Skip to content

Commit 23a5119

Browse files
ppham27sayakpaul
authored andcommitted
Check shape and remove deprecated APIs in scheduling_ddpm_flax.py (#7703)
`model_output.shape` may only have rank 1. There are warnings related to use of random keys. ``` tests/schedulers/test_scheduler_flax.py: 13 warnings /Users/phillypham/diffusers/src/diffusers/schedulers/scheduling_ddpm_flax.py:268: FutureWarning: normal accepts a single key, but was given a key array of shape (1, 2) != (). Use jax.vmap for batching. In a future JAX version, this will be an error. noise = jax.random.normal(split_key, shape=model_output.shape, dtype=self.dtype) tests/schedulers/test_scheduler_flax.py::FlaxDDPMSchedulerTest::test_betas /Users/phillypham/virtualenv/diffusers/lib/python3.9/site-packages/jax/_src/random.py:731: FutureWarning: uniform accepts a single key, but was given a key array of shape (1,) != (). Use jax.vmap for batching. In a future JAX version, this will be an error. u = uniform(key, shape, dtype, lo, hi) # type: ignore[arg-type] ```
1 parent c8a4893 commit 23a5119

File tree

1 file changed

+7
-3
lines changed

1 file changed

+7
-3
lines changed

src/diffusers/schedulers/scheduling_ddpm_flax.py

Lines changed: 7 additions & 3 deletions
Original file line numberDiff line numberDiff line change
@@ -222,9 +222,13 @@ def step(
222222
t = timestep
223223

224224
if key is None:
225-
key = jax.random.PRNGKey(0)
225+
key = jax.random.key(0)
226226

227-
if model_output.shape[1] == sample.shape[1] * 2 and self.config.variance_type in ["learned", "learned_range"]:
227+
if (
228+
len(model_output.shape) > 1
229+
and model_output.shape[1] == sample.shape[1] * 2
230+
and self.config.variance_type in ["learned", "learned_range"]
231+
):
228232
model_output, predicted_variance = jnp.split(model_output, sample.shape[1], axis=1)
229233
else:
230234
predicted_variance = None
@@ -264,7 +268,7 @@ def step(
264268

265269
# 6. Add noise
266270
def random_variance():
267-
split_key = jax.random.split(key, num=1)
271+
split_key = jax.random.split(key, num=1)[0]
268272
noise = jax.random.normal(split_key, shape=model_output.shape, dtype=self.dtype)
269273
return (self._get_variance(state, t, predicted_variance=predicted_variance) ** 0.5) * noise
270274

0 commit comments

Comments
 (0)