
    uki(                        d dl mZmZ d dlZd dlmZ d dlZd dlm	Z	 d dlm
Z
 d dlmZ d dlmZ d dlmZ d d	lmZ d d
lmZ d dlmZ d dlmZ d dlmZ  ej*                  d      Z eej.                        Z eej0                        Z eej2                        Zdddddded   dedededz  deeef   dz  defdZ e	j@                  d      Z!de!_"         e
jF                  e!        G d dejH                        Z% e%       Z&ejN                  jQ                  e%       ejR                  jQ                  e%       e!jT                  de+e	jX                  df   defd       Z-d Z.e.ej^                  e!<   d  Z0e0ejb                  e!<    ejd                  ejf                  e!      ejh                  e!<   d!ejj                  d"ed#ed$ededee+eef      defd%Z6 ejn                  e!e6       y)&    )CallableSequenceN)Any)core)dispatch)effects)ffi)	tree_util)util)ad)batching)mlirz jax.experimental.buffer_callbackF)has_side_effectvmap_methodinput_output_aliasescommand_buffer_compatiblecallback).Nresult_shape_dtypesr   r   r   r   c                ~    	 t        j                  |      \  }	t        d |D               	fd}|S )a  An experimental callback that operates in place on device buffers.

  Only supported on CPU and GPU backends.

  Note that the plan is for this to eventually be replaced by a consolidated
  callback API built using JAX mutable arrays, but for now this provides a
  mechanism for prototyping computational kernels using other Python libraries
  including Numpy, PyTorch, Cupy, and others.

  Let's start with a simple example:

    >>> def py_add_one_inplace(ctx, out, x):
    ...   np.asarray(out)[...] = np.asarray(x) + 1
    ...
    >>> x = jnp.array(41, dtype=jnp.int32)
    >>> out_type = jax.ShapeDtypeStruct(x.shape, x.dtype)
    >>> add_one = buffer_callback(py_add_one_inplace, out_type)
    >>> add_one(x)  # doctest: +SKIP
    Array(42, dtype=int32)

  In this example, we're executing a numpy computation via JAX, and this could
  have been implemented using :func:`jax.pure_callback`, but in this case, the
  output is being populated in-place. This means that JAX doesn't need to copy
  the output arrays upon returning from the callback. Note that even though the
  callback function operates on mutable buffers, JAX still sees this as an
  operation that consumes and produces regular immutable JAX arrays.

  Unlike the other JAX callback APIs, ``buffer_callback`` requires that the
  user-defined Python function have the following signature:

  .. code-block:: python

    def callback(ctx: ExecutionContext, out, *args) -> None:
      ...

  where ``ctx`` is an instance of
  :class:`~jax.experimental.buffer_callback.ExecutionContext`, which mainly
  provides access to XLA's computation stream when running on GPU, ``out`` is a
  pytree of mutable :class:`~jax.experimental.buffer_callback.Buffer` objects,
  and the ``args`` arguments have the same pytree structure as the inputs, but
  each leaf is :class:`~jax.experimental.buffer_callback.Buffer`. This callback
  should not return any values, and it should overwrite the ``out`` buffers in
  place to output values back to JAX.

  It's important to note that this Python function can't really be called
  except via ```buffer_callback`` itself, because it's not (yet!) possible to
  construct mutable JAX buffers directly in Python.

  The bespoke :class:`~jax.experimental.buffer_callback.Buffer` type is an
  array-like object that supports the ``__array__`` protocol on CPU, the
  ``__cuda_array_interface__`` protocol on GPU, and the ``__dlpack__`` protocol
  on both CPU and GPU.

  Args:
    callback: A Python function with the signature and behavior described above.
    result_shape_dtypes: A pytree whose leaves have ``shape`` and ``dtype``
      attributes, with a structure that matches the expected output of the
      callback function at runtime. :class:`jax.ShapeDtypeStruct` is often used
      to define leaf values.
    has_side_effect: Whether the callback has side effects.
    vmap_method: A string specifying how the callback transforms under
      :func:`~jax.vmap` as described in the docs for :func:`~jax.pure_callback`.
    input_output_aliases: a dictionary mapping the index of some inputs to
      the index of the output that aliases them. These indices are in the
      flattened inputs and outputs.
    command_buffer_compatible: if ``True``, the callback will be traced into
      the command buffer. This means that the Python code should only be
      executed once, and then the operations will be replayed for every
      subsequent call.

  Returns:
    A new callable that accepts :class:`jax.Array` inputs (and pytrees thereof),
    and  pytree of :class:`jax.Array` objects whose structure matches that
    of ``result_shape_dtypes``.

  See Also:
    - :func:`jax.pure_callback`: callback designed for pure host functions.
    - :func:`jax.experimental.io_callback`: callback designed for impure host
      functions.
    - :func:`jax.debug.callback`: callback designed for general-purpose
      debugging.
    - :func:`jax.debug.print`: callback designed for printing.
  c              3   p   K   | ].  }t        j                  |j                  |j                         0 y wN)r   ShapedArrayshapedtype).0xs     S/home/cdr/jupyterlab/.venv/lib/python3.12/site-packages/jax/_src/buffer_callback.py	<genexpr>z"buffer_callback.<locals>.<genexpr>   s*      -.dqww(s   46c                     t        j                  | |f      \  }}|D cg c]  }t        j                  |       }}d}t	        j                               D ]  \  }}t        |      t        |      }}|t        |       k\  r!t        d| d| d| dt        |        d	      |t              k\  r!t        d| d| d| dt               d	      ||   }	|   }
t        j                  |	|
      st        d| d| d|	 d	|
 d
	      |||ffz  } t        j                  |||d}t        j                  |      S c c}w )N z+input_output_aliases contains the mapping ':z' with input index z outside the range [0, z).z' with output index z,' referring to an input with abstract value z/ and an output with a different abstract value .)r   result_avalsin_treeout_treer   r   r   r   )r
   tree_flattenr   get_avalsorteditemsintlen
ValueErrorr	   _check_compatible_avalsbuffer_callback_pbindtree_unflatten)argskwargs	flat_argsr$   r   in_avalsstatic_input_output_aliasesi_idxo_idxin_avalout_avalout_flatr   r   flat_result_avalsr   r   r%   r   s               r   wrapped_callbackz)buffer_callback.<locals>.wrapped_callback   s   "//v?Iw*34Qa 4H4?A' !5!;!;!=> 9,%5z3u:uCI;E7!E7 K""'(?T2    C)**;E7!E7 K##(')@&'(,- - 5/$U+**7H=;E7!E7 K;;B) D88@zDE E 	$'88#'9* !%%	&'8";
H ##Hh77G 5s   E	)r
   r&   tuple)
r   r   r   r   r   r   flat_shape_dtypesr<   r;   r%   s
   ` ````  @@r   buffer_callbackr?   &   sH    x !* 6 67J KX 2C &8 &8P 
    r?   Tc                       e Zd Zd Zy)BufferCallbackEffectc                      y)NBufferCallbackr    )selfs    r   __str__zBufferCallbackEffect.__str__   s    r@   N)__name__
__module____qualname__rF   r    r@   r   rB   rB      s    r@   rB   r#   .c                 >    ~|rt         hnt        j                  }| |fS r   )_BufferCallbackEffectr   
no_effects)r#   r   r1   _r   s        r   _buffer_callback_abstract_evalrN      s#     '6"#DOO'	w	r@   c                      ~ ~t        d      )NziBuffer callbacks do not support JVP. Please use `jax.custom_jvp` to use callbacks while taking gradients.r,   r1   r2   s     r   _buffer_callback_jvp_rulerR          
FM	N Nr@   c                      ~ ~t        d      )NzoBuffer callbacks do not support transpose. Please use `jax.custom_vjp` to use callbacks while taking gradients.rP   rQ   s     r   _buffer_callback_transpose_rulerU      rS   r@   ctxr1   r$   r%   c                   t        | j                  j                        dkD  rt        d      | j                  j                  d   }	ddddj	                  |	      }
|
t        d|	 d      |r	|	d	v r|
d
z  }
dt        ffd}| j                  j                  |       t        j                  t        | j                  j                        dz
        }t        j                  |
|t        |            } || g|d|iS )N   z+multi-platform lowering for buffer_callbackr   xla_buffer_python_cpu_callbackxla_buffer_python_gpu_callback)cpucudarocmz#`buffer_callback` not supported on z	 backend.)r\   r]   _cmd_bufferr1   c                     t        j                  |j                  g      \  }}t        j                  |      \  }}t        j                  	|      } | |g|i |t        d      y)Nz4buffer_callback callback must not return any values.r    )r   
split_list
num_leavesr
   r0   r,   )
exec_ctxr1   args_inargs_out
py_args_inpy_kwargs_inpy_args_outr   r$   r%   s
          r   r<   z3_buffer_callback_lowering.<locals>.wrapped_callback   sp    w/A/A.BCGX(77IJ**8X>K+C
ClCOMNNr@   )r   operand_output_aliasesindex)r+   module_context	platformsNotImplementedErrorgetr,   r   add_host_callbacknpuint64host_callbacksr	   ffi_loweringdict)rV   r   r$   r%   r   r   r   r1   rM   platformtarget_namer<   ri   rules    ```          r   _buffer_callback_loweringrw      s    				%	%&*
K
LL))!,(-.. CM	 
 
:8*IN
OO8/?#?= K  &&'78
))C**99:Q>
?%			%!"67
$
 
c	&D	&	&&r@   )8collections.abcr   r   	functoolstypingr   numpyro   jax._srcr   r   r   r	   r
   r   jax._src.interpretersr   r   r   jax._src.libffi_lib
set_moduleexportBufferExecutionStageExecutionContextobjectboolstrrs   r*   r?   	Primitiver.   multiple_resultssimple_implEffectrB   rK   lowerable_effectsadd_typecontrol_flow_allowed_effectsdef_effectful_abstract_evalr=   r   rN   rR   primitive_jvpsrU   primitive_transposespartialffi_batching_ruleprimitive_batchersLoweringRuleContextrw   register_loweringr    r@   r   <module>r      sX   /          $ * & '	;	<		../'223  ""26&+Iy!II 	I
 tI sCx.4/I  $IX #DNN#45 %)  "   & '7>>  -.     " "#7 8  $ $ - -.B C ..((#-.  /N
 (A  # $N
 .M  ) *1B1B1B,2  - .
)'		!	!)')' 	)'
 )' )' #5c?3)'  $)'T   (*C Dr@   