Skip to content

Commit

Permalink
feat(training-utils): support device and dtype params in compute_dens…
Browse files Browse the repository at this point in the history
…ity_for_timestep_sampling (#10699)

* feat(training-utils): support device and dtype params in compute_density_for_timestep_sampling

* chore: update type hint

* refactor: use union for type hint

---------

Co-authored-by: Sayak Paul <[email protected]>
  • Loading branch information
badayvedat and sayakpaul authored Feb 1, 2025
1 parent 5d2d239 commit 9f28f1a
Showing 1 changed file with 10 additions and 5 deletions.
15 changes: 10 additions & 5 deletions src/diffusers/training_utils.py
Original file line number Diff line number Diff line change
Expand Up @@ -248,7 +248,13 @@ def _set_state_dict_into_text_encoder(


def compute_density_for_timestep_sampling(
weighting_scheme: str, batch_size: int, logit_mean: float = None, logit_std: float = None, mode_scale: float = None
weighting_scheme: str,
batch_size: int,
logit_mean: float = None,
logit_std: float = None,
mode_scale: float = None,
device: Union[torch.device, str] = "cpu",
generator: Optional[torch.Generator] = None,
):
"""
Compute the density for sampling the timesteps when doing SD3 training.
Expand All @@ -258,14 +264,13 @@ def compute_density_for_timestep_sampling(
SD3 paper reference: https://arxiv.org/abs/2403.03206v1.
"""
if weighting_scheme == "logit_normal":
# See 3.1 in the SD3 paper ($rf/lognorm(0.00,1.00)$).
u = torch.normal(mean=logit_mean, std=logit_std, size=(batch_size,), device="cpu")
u = torch.normal(mean=logit_mean, std=logit_std, size=(batch_size,), device=device, generator=generator)
u = torch.nn.functional.sigmoid(u)
elif weighting_scheme == "mode":
u = torch.rand(size=(batch_size,), device="cpu")
u = torch.rand(size=(batch_size,), device=device, generator=generator)
u = 1 - u - mode_scale * (torch.cos(math.pi * u / 2) ** 2 - 1 + u)
else:
u = torch.rand(size=(batch_size,), device="cpu")
u = torch.rand(size=(batch_size,), device=device, generator=generator)
return u


Expand Down

0 comments on commit 9f28f1a

Please sign in to comment.