jax.lax.ge
Contents
jax.lax.ge#
- jax.lax.ge(x,y)[source]#
Elementwise greater-than-or-equals:\(x \geq y\).
This function lowers directly to thestablehlo.compare operationwith
comparison_direction=GEandcompare_typeset accordingto the input dtype.- Parameters:
x (ArrayLike) – Input arrays. Must have matching non-complex dtypes. If neither isa scalar,
xandymust have the same number of dimensions andbe broadcast compatible.y (ArrayLike) – Input arrays. Must have matching non-complex dtypes. If neither isa scalar,
xandymust have the same number of dimensions andbe broadcast compatible.
- Returns:
A boolean array of shape
lax.broadcast_shapes(x.shape,y.shape)containing the elementwise greater-than-or-equal comparison.- Return type:
See also
jax.numpy.greater_equal(): NumPy wrapper for this API, alsoaccessible via thex>=yoperator on JAX arrays.jax.lax.eq(): elementwise equaljax.lax.ne(): elementwise not-equaljax.lax.gt(): elementwise greater-thanjax.lax.le(): elementwise less-than-or-equaljax.lax.lt(): elementwise less-than
