Skip to content

misc fixes #2282

New issue

Have a question about this project? Sign up for a free GitHub account to open an issue and contact its maintainers and the community.

By clicking “Sign up for GitHub”, you agree to our terms of service and privacy statement. We’ll occasionally send you account related emails.

Already on GitHub? Sign in to your account

Merged
merged 2 commits into from
Feb 8, 2023
Merged
Show file tree
Hide file tree
Changes from all commits
Commits
File filter

Filter by extension

Filter by extension

Conversations
Failed to load comments.
Loading
Jump to
Jump to file
Failed to load files.
Loading
Diff view
Diff view
11 changes: 7 additions & 4 deletions scripts/convert_original_stable_diffusion_to_diffusers.py
Original file line number Diff line number Diff line change
Expand Up @@ -48,7 +48,10 @@
"--pipeline_type",
default=None,
type=str,
help="The pipeline type. If `None` pipeline will be automatically inferred.",
help=(
"The pipeline type. One of 'FrozenOpenCLIPEmbedder', 'FrozenCLIPEmbedder', 'PaintByExample'"
". If `None` pipeline will be automatically inferred."
),
)
parser.add_argument(
"--image_size",
Expand All @@ -65,7 +68,7 @@
type=str,
help=(
"The prediction type that the model was trained on. Use 'epsilon' for Stable Diffusion v1.X and Stable"
" Siffusion v2 Base. Use 'v-prediction' for Stable Diffusion v2."
" Diffusion v2 Base. Use 'v_prediction' for Stable Diffusion v2."
Copy link
Contributor Author

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

elif self.config.prediction_type == "v_prediction":

The scheduler takes 'v_prediction'

Copy link
Contributor

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Good catch!

),
)
parser.add_argument(
Expand All @@ -79,8 +82,7 @@
)
parser.add_argument(
"--upcast_attention",
default=False,
type=bool,
action="store_true",
help=(
"Whether the attention computation should always be upcasted. This is necessary when running stable"
" diffusion 2.1."
Expand Down Expand Up @@ -111,5 +113,6 @@
num_in_channels=args.num_in_channels,
upcast_attention=args.upcast_attention,
from_safetensors=args.from_safetensors,
device=args.device,
)
pipe.save_pretrained(args.dump_path, safe_serialization=args.to_safetensors)
7 changes: 1 addition & 6 deletions src/diffusers/pipelines/audio_diffusion/mel.py
Original file line number Diff line number Diff line change
Expand Up @@ -13,17 +13,12 @@
# limitations under the License.


import warnings
import numpy as np # noqa: E402

from ...configuration_utils import ConfigMixin, register_to_config
from ...schedulers.scheduling_utils import SchedulerMixin


warnings.filterwarnings("ignore")
Copy link
Contributor Author

@williamberman williamberman Feb 8, 2023

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

This warning filter will apply globally because this module ends up being imported at diffusers/__init__.py


import numpy as np # noqa: E402


try:
import librosa # noqa: E402

Expand Down
12 changes: 9 additions & 3 deletions src/diffusers/pipelines/stable_diffusion/convert_from_ckpt.py
Original file line number Diff line number Diff line change
Expand Up @@ -39,10 +39,13 @@
from diffusers.pipelines.paint_by_example import PaintByExampleImageEncoder, PaintByExamplePipeline
from diffusers.pipelines.stable_diffusion import StableDiffusionSafetyChecker

from ...utils import is_omegaconf_available, is_safetensors_available
from ...utils import is_omegaconf_available, is_safetensors_available, logging
from ...utils.import_utils import BACKENDS_MAPPING


logger = logging.get_logger(__name__) # pylint: disable=invalid-name


def shave_segments(path, n_shave_prefix_segments=1):
"""
Removes segments. Positive values shave the first segments, negative shave the last segments.
Expand Down Expand Up @@ -801,11 +804,11 @@ def load_pipeline_from_original_stable_diffusion_ckpt(
corresponding to the original architecture. If `None`, will be
automatically inferred by looking for a key that only exists in SD2.0 models.
:param image_size: The image size that the model was trained on. Use 512 for Stable Diffusion v1.X and Stable
Siffusion v2
Diffusion v2
Base. Use 768 for Stable Diffusion v2.
:param prediction_type: The prediction type that the model was trained on. Use `'epsilon'` for Stable Diffusion
v1.X and Stable
Siffusion v2 Base. Use `'v-prediction'` for Stable Diffusion v2.
Diffusion v2 Base. Use `'v_prediction'` for Stable Diffusion v2.
:param num_in_channels: The number of input channels. If `None` number of input channels will be automatically
inferred. :param scheduler_type: Type of scheduler to use. Should be one of `["pndm", "lms", "heun", "euler",
"euler-ancestral", "dpm", "ddim"]`. :param model_type: The pipeline type. `None` to automatically infer, or one of
Expand All @@ -820,6 +823,8 @@ def load_pipeline_from_original_stable_diffusion_ckpt(
`checkpoint_path` is in `safetensors` format, load checkpoint with safetensors instead of PyTorch. :return: A
StableDiffusionPipeline object representing the passed-in `.ckpt`/`.safetensors` file.
"""
if prediction_type == "v-prediction":
prediction_type = "v_prediction"
Comment on lines +826 to +827
Copy link
Contributor Author

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Just in case because of previous script docs


if not is_omegaconf_available():
raise ValueError(BACKENDS_MAPPING["omegaconf"][1])
Expand Down Expand Up @@ -957,6 +962,7 @@ def load_pipeline_from_original_stable_diffusion_ckpt(
# Convert the text model.
if model_type is None:
model_type = original_config.model.params.cond_stage_config.target.split(".")[-1]
logger.debug(f"no `model_type` given, `model_type` inferred as: {model_type}")

if model_type == "FrozenOpenCLIPEmbedder":
text_model = convert_open_clip_checkpoint(checkpoint)
Expand Down
1 change: 0 additions & 1 deletion src/diffusers/schedulers/scheduling_ddim.py
Original file line number Diff line number Diff line change
Expand Up @@ -305,7 +305,6 @@ def step(
prev_sample = alpha_prod_t_prev ** (0.5) * pred_original_sample + pred_sample_direction

if eta > 0:
# randn_like does not support generator https://github.com/pytorch/pytorch/issues/27072
Copy link
Contributor Author

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

left over comment before we switched to rand_tensor

device = model_output.device
if variance_noise is not None and generator is not None:
raise ValueError(
Expand Down
5 changes: 4 additions & 1 deletion src/diffusers/schedulers/scheduling_unclip.py
Original file line number Diff line number Diff line change
Expand Up @@ -106,8 +106,11 @@ def __init__(
clip_sample: bool = True,
clip_sample_range: Optional[float] = 1.0,
prediction_type: str = "epsilon",
beta_schedule: str = "squaredcos_cap_v2",
):
# beta scheduler is "squaredcos_cap_v2"
if beta_schedule != "squaredcos_cap_v2":
raise ValueError("UnCLIPScheduler only supports `beta_schedule`: 'squaredcos_cap_v2'")

self.betas = betas_for_alpha_bar(num_train_timesteps)

self.alphas = 1.0 - self.betas
Expand Down
29 changes: 21 additions & 8 deletions src/diffusers/utils/testing_utils.py
Original file line number Diff line number Diff line change
Expand Up @@ -17,23 +17,36 @@
from packaging import version

from .import_utils import is_flax_available, is_onnx_available, is_torch_available
from .logging import get_logger


global_rng = random.Random()

logger = get_logger(__name__)

if is_torch_available():
import torch

torch_device = "cuda" if torch.cuda.is_available() else "cpu"
is_torch_higher_equal_than_1_12 = version.parse(version.parse(torch.__version__).base_version) >= version.parse(
"1.12"
)
if "DIFFUSERS_TEST_DEVICE" in os.environ:
torch_device = os.environ["DIFFUSERS_TEST_DEVICE"]

if is_torch_higher_equal_than_1_12:
# Some builds of torch 1.12 don't have the mps backend registered. See #892 for more details
mps_backend_registered = hasattr(torch.backends, "mps")
torch_device = "mps" if (mps_backend_registered and torch.backends.mps.is_available()) else torch_device
available_backends = ["cuda", "cpu", "mps"]
if torch_device not in available_backends:
raise ValueError(
f"unknown torch backend for diffusers tests: {torch_device}. Available backends are:"
f" {available_backends}"
)
logger.info(f"torch_device overrode to {torch_device}")
Comment on lines +30 to +39
Copy link
Contributor Author

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

It can be helpful to force the tests to run on a particular device -- i.e. I use this snippet when I want to force the tests to run on the cpu when I'm using a macbook or a machine with cuda

Copy link
Contributor

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Looks good to me!

else:
torch_device = "cuda" if torch.cuda.is_available() else "cpu"
is_torch_higher_equal_than_1_12 = version.parse(
version.parse(torch.__version__).base_version
) >= version.parse("1.12")

if is_torch_higher_equal_than_1_12:
# Some builds of torch 1.12 don't have the mps backend registered. See #892 for more details
mps_backend_registered = hasattr(torch.backends, "mps")
torch_device = "mps" if (mps_backend_registered and torch.backends.mps.is_available()) else torch_device


def torch_all_close(a, b, *args, **kwargs):
Expand Down
1 change: 1 addition & 0 deletions tests/pipelines/unclip/test_unclip.py
Original file line number Diff line number Diff line change
Expand Up @@ -30,6 +30,7 @@

class UnCLIPPipelineFastTests(PipelineTesterMixin, unittest.TestCase):
pipeline_class = UnCLIPPipeline
test_xformers_attention = False

required_optional_params = [
"generator",
Expand Down
1 change: 1 addition & 0 deletions tests/test_pipelines_common.py
Original file line number Diff line number Diff line change
Expand Up @@ -259,6 +259,7 @@ def _test_inference_batch_single_identical(
# Taking the median of the largest <n> differences
# is resilient to outliers
diff = np.abs(output_batch[0][0] - output[0][0])
diff = diff.flatten()
diff.sort()
max_diff = np.median(diff[-5:])
else:
Expand Down