Source code for blackjax.util
"""Utility functions for BlackJax."""
from functools import partial
from typing import Union
import jax.numpy as jnp
from jax import jit, lax
from jax.flatten_util import ravel_pytree
from jax.random import normal
from jax.tree_util import tree_leaves
from blackjax.types import Array, ArrayLikeTree, ArrayTree, PRNGKey
@partial(jit, static_argnames=("precision",), inline=True)
[docs]def linear_map(diag_or_dense_a, b, *, precision="highest"):
"""Perform a linear map of the form y = Ax.
Dispatch matrix multiplication to either jnp.dot or jnp.multiply.
Unlike jax.numpy.dot, this function output an Array that match the dtype
and shape of the 2nd input:
- diag_or_dense_a is a scalar or 1d vector, `diag_or_dense_a * b` is returned
- diag_or_dense_a is a 2d matrix, `diag_or_dense_a @ b` is returned
Note that unlike jax.numpy.dot, here we defaults to full (highest)
precision. This is more useful for numerical algorithms and will be the
default for jax.numpy in the future:
https://github.com/google/jax/pull/7859
Parameters
----------
diag_or_dense_a:
A diagonal (1d vector) or dense matrix (2d square matrix).
b:
A vector.
precision:
The precision of the computation. See jax.lax.dot_general for
more details.
Returns
-------
The result vector of the matrix multiplication.
"""
dtype = jnp.result_type(diag_or_dense_a.dtype, b.dtype)
diag_or_dense_a = diag_or_dense_a.astype(dtype)
b = b.astype(dtype)
ndim = jnp.ndim(diag_or_dense_a)
if ndim <= 1:
return lax.mul(diag_or_dense_a, b)
else:
return lax.dot(diag_or_dense_a, b, precision=precision)
# TODO(https://github.com/blackjax-devs/blackjax/issues/376)
# Refactor this function to not use ravel_pytree might be more performant.
[docs]def generate_gaussian_noise(
rng_key: PRNGKey,
position: ArrayLikeTree,
mu: Union[float, Array] = 0.0,
sigma: Union[float, Array] = 1.0,
) -> ArrayTree:
"""Generate N(mu, sigma) noise with output structure that match a given PyTree.
Parameters
----------
rng_key:
The pseudo-random number generator key used to generate random numbers.
position:
PyTree that the structure the output should to match.
mu:
The mean of the Gaussian distribution.
sigma:
The standard deviation of the Gaussian distribution.
Returns
-------
Gaussian noise following N(mu, sigma) that match the structure of position.
"""
p, unravel_fn = ravel_pytree(position)
sample = normal(rng_key, shape=p.shape, dtype=p.dtype)
return unravel_fn(mu + linear_map(sigma, sample))
[docs]def pytree_size(pytree: ArrayLikeTree) -> int:
"""Return the dimension of the flatten PyTree."""
return sum(jnp.size(value) for value in tree_leaves(pytree))
[docs]def index_pytree(input_pytree: ArrayLikeTree) -> ArrayTree:
"""Builds a PyTree with elements indicating its corresponding index on a flat array.
Various algorithms in BlackJAX take as input a 1 or 2 dimensional array which somehow
affects the sampling or approximation of a PyTree. For instance, in HMC a 1 or 2
dimensional inverse mass matrix is used when simulating Hamilonian dynamics on
PyTree position and momentum variables. It is usually unclear how the elements of the
array interact with the PyTree. This function demonstrates how all algorithms map an
array to a PyTree of equivalent dimension.
The function returns the index of a 1 dimensional array corresponding to each element of
the PyTree. This way the user can tell which element in the PyTree corresponds to which
column (and row) of a 1 dimensional (or 2 dimensional) array.
Parameters
----------
input_pytree:
Example PyTree.
Returns
-------
PyTree mapping each individual element of an arange array to elements in the PyTree.
"""
flat_input, unravel_fn = ravel_pytree(input_pytree)
(dim_input,) = flat_input.shape
array = jnp.arange(dim_input, dtype=flat_input.dtype)
return unravel_fn(array)