Skip to content

Commit

Permalink
Add ChiSquareRV JAX implementation
Browse files Browse the repository at this point in the history
  • Loading branch information
larryshamalama authored and brandonwillard committed Mar 9, 2023
1 parent 83ce997 commit 2fa3f58
Show file tree
Hide file tree
Showing 2 changed files with 28 additions and 1 deletion.
16 changes: 15 additions & 1 deletion aesara/link/jax/dispatch/random.py
Original file line number Diff line number Diff line change
Expand Up @@ -370,7 +370,6 @@ def jax_sample_fn_wald(op):
def sample_fn(rng, size, dtype, *parameters):
rng_key = rng["jax_state"]
rng_key, sampling_key = jax.random.split(rng_key, 2)

mean, scale = parameters

key1, key2 = jax.random.split(sampling_key, 2)
Expand All @@ -391,6 +390,21 @@ def sample_fn(rng, size, dtype, *parameters):
return sample_fn


@jax_sample_fn.register(aer.ChiSquareRV)
def jax_sample_fn_chisquare(op):
"""JAX implementation of `ChiSquareRV`"""

def sample_fn(rng, size, dtype, *parameters):
rng_key = rng["jax_state"]
rng_key, sampling_key = jax.random.split(rng_key, 2)
(df,) = parameters
sample = jax.random.gamma(sampling_key, df / 2, size, dtype) * 2
rng["jax_state"] = rng_key
return (rng, sample)

return sample_fn


@jax_sample_fn.register(aer.GeometricRV)
def jax_sample_fn_geometric(op):
"""JAX implementation of `GeometricRV`."""
Expand Down
13 changes: 13 additions & 0 deletions tests/link/jax/test_random.py
Original file line number Diff line number Diff line change
Expand Up @@ -97,6 +97,19 @@ def test_random_updates(rng_ctor):
lambda *args: args,
None,
),
(
aer.chisquare,
[
set_test_value(
at.dvector(),
np.array([1.0, 2.0], dtype=np.float64),
)
],
(2,),
"chi2",
lambda *args: args,
None,
),
(
aer.exponential,
[
Expand Down

0 comments on commit 2fa3f58

Please sign in to comment.