
    uki>D                        d dl mZ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 d dlmZ d dlmZ d dlmZ d dlmZ  ej.                  e        ej.                  e       ej2                  ecZZej8                  ecZZe	j>                   G d d             Z d Z!ejD                  dedef   dejF                  de$de$dee%   deejL                     fd       Z'dejP                  fdZ)dejP                  fdZ*dejV                  d e,dejP                  d!ede-ejP                  ee%   f   f   fd"Z.dejP                  fd#Z/dee%   d$ej`                  fd%Z1 ejd                  d&      Z3d'e3_4        e3jk                  e)       e3jm                  e*        ejn                  e3 ejp                  e)d'(             e.ejr                  e3<   e/ejt                  e3<   e1ejv                  e3<   d)ejx                  d*ejd                  d+ejd                  d,ejx                  fd-Z= ejd                  d.      Z>d'e>_4        e>jk                  e)       e>jm                  e*       y)/    )CallableSequenceN)Any)api_util)core)custom_api_util)errors)linear_util)source_info_util)traceback_util)	tree_util)util)ad)batching)mlir)partial_evalc                       e Zd ZU dZedef   ed<   ee   ed<   edef   dz  ed<   dddedef   dee   fd	Z	e
j                  Zdedef   d
edef   fdZej                  d        Zy)
custom_dcea  Customize the DCE behavior of a JAX-transformable function.

  JAX uses dead code elimination (DCE) to remove unused computations from a
  JAX program. This typically works transparently when the program is
  completely specified by known JAX operations, but opaque kernels like calls
  to :py:func:`~jax.experimental.pallas.pallas_call` or
  :py:func:`~jax.ffi.ffi_call`, for example, may cause problems.

  In JAX, DCE is performed when a function is staged out using
  :py:func:`jax.jit`, so it won't be applied when running JAX in eager mode.
  Similarly, the ``custom_dce`` decorator requires that both the decorated
  function and the custom DCE rule be compatible with :py:func:`~jax.jit`.

  This decorator allows users to customize the DCE behavior of a function by
  defining a custom DCE rule. For a ``custom_dce`` wrapped function
  ``f(*args)``, the signature of the DCE rule is ``dce_rule(used_outs, *args)``
  where ``used_outs`` is a Pytree with the same structure as the output of
  ``f``, and each leaf is is a ``bool`` indicating which outputs should
  be computed. The remaining arguments ``*args`` are the original arguments to
  ``f``. The rule ``dce_rule`` should return a Pytree with the same structure
  as the original output of ``f``, but any unused outputs can be replaced with
  ``None``.

  For example::

    >>> @jax.experimental.custom_dce.custom_dce
    ... def f(x, y):
    ...   return jnp.sin(x) * y, x * jnp.sin(y)
    ...
    >>> @f.def_dce
    ... def f_dce_rule(used_outs, x, y):
    ...   return (
    ...       jnp.sin(x) * y if used_outs[0] else None,
    ...       x * jnp.sin(y) if used_outs[1] else None,
    ...   )

  In this example, ``used_outs`` is a ``tuple`` with two ``bool`` values,
  indicating which outputs are required. The DCE rule only computes the
  required outputs, replacing the unused outputs with ``None``.

  If the ``static_argnums`` argument is provided to ``custom_dce``, the
  indicated arguments are treated as static when the function is traced, and
  they will be moved to the front when calling the DCE rule. For example, if
  ``fun`` takes 2 arguments ``fun(x, y)``, and ``static_argnums`` is ``(1,)``,
  then the DCE rule will be called as ``dce_rule(y, used_outs, x)``.
  .funstatic_argnumsNdce_rule r   c                Z    t        j                  | |       || _        || _        d | _        y N)	functoolsupdate_wrapperr   r   r   )selfr   r   s      N/home/cdr/jupyterlab/.venv/lib/python3.12/site-packages/jax/_src/custom_dce.py__init__zcustom_dce.__init__]   s*     T3'DH(DDM    returnc                     || _         |S )a2  Define a custom DCE rule for this function.

    Args:
      dce_rule: A function that takes (a) any arguments indicated as static
        using ``static_argnums``, (b) a Pytree of ``bool`` values
        (``used_outs``) indicating which outputs should be computed, and (c)
        the rest of the (non-static) arguments to the original function. The
        rule should return a Pytree with with the same structure as the output
        of the original function, but any unused outputs (as indicated by
        ``used_outs``) can be replaced with ``None``.
    )r   )r   r   s     r   def_dcezcustom_dce.def_dceg   s     DMOr!   c           
          t        j                   j                         j                  t	        d d      t        j                   j                        t        j                  d j                  |i  j                        }t        j                  d j                  |i  j                        }	 t        j                   j                  ||      } j                  rt         j                        }|D ]  }t        ||           t        t        |            D cg c]	  }||vs| }}t        j                  t!        j"                   j                  |      ||d	
      \  }	}
 j                  D cg c]  }||   	 }}t        j$                  t!        j"                   j                  |      |      nDt!        j"                   j                  |      }	t!        j"                   j                  |      |}
t'        j(                  |
      \  }t        j*                  |	      \  }|D cg c]  }t-        j.                  |       c}t0        j2                  dt4        dt6        t,        j8                  t:        t4           f   f fd       }t1        j<                  |      \  }}}t1        j>                  t1        j@                  |            }|jB                  tE        jF                  g ||t        |      ||d}t'        jH                          |      S # t        $ r!}t        d|j                   d|       |d }~ww xY wc c}w c c}w c c}w )Nz,No DCE rule defined for custom_dce function z using def_dce.r   r   custom_dce_rulez9The input arguments to the custom_dce-decorated function zT could not be resolved to positional-only arguments. Binding failed with the error:

debug_infoF)require_static_args_hashable	used_outsr"   c            	         j                   D ]  }|s|j                           t        	|               \  }}j                  J t	        j
                  |
      \  }}}t        |t        t              }t	        j                  |dgt        |j                        z        \  }}t        |t        t              }t        j                  ||      |fS )NT)storesresetflatten_dce_ruler   petrace_to_jaxpr_dynamicswap_primitivescustom_dce_pdce_sential_p	dce_jaxprlenoutvarsr   ClosedJaxpr)r*   store	flat_rulerule_out_treer4   _
dce_constsused_insr   fun_namein_avalsin_tree	out_avalsout_tree	rule_namer   s           r   dce_jaxpr_thunkz,custom_dce.__call__.<locals>.dce_jaxpr_thunk   s     ?? %
++- "2





*
"i ]]&&&!#!:!:
X"iJ ")\=IiLL
dVc)"3"344i ")]LIii4h>>r!   
num_consts	fun_jaxprrD   )%r   r>   r   r   AttributeErrorr   r(   r   resolve_kwargs	TypeError	func_namesetcheck_for_tracersranger5   argnums_partiallu	wrap_initprepend_static_argsr   tree_flattenflatten_fun_nokwargsr   get_avalr/   _memoizebooltupler7   r   r0   close_jaxprconvert_constvars_jaxprrA   r2   bindtree_unflatten)r   argskwargsdebug
debug_ruleer   idyn_argnumsr   dyn_argsstatic_args	args_flatflat_funxrD   jaxprr;   constsclosed_callout_flatr   r>   r?   r@   rA   rB   rC   s   `                    @@@@@@@r   __call__zcustom_dce.__call__y   s   }}TXX&H}}8
 C   dmm,Idhh $b/3/B/BDE $$%6%)2484G4GIJ$$TXXtV<d 4../n #!$q'"# %c$i 0L1A^4KQLkL..
,,txxE
2

',	mc8 '+&9&9:T!W:k:--
,,t}}
<kh LLe4cdmm
Chh"//9Iw!66sGDHh*34Qa 4H[[??	t$/	0? ? ?@ 008DE1f..!;!;E!BCK%%I   		 v;'H ##HJ99U  
E__ 889s< 	 M ; 5s0   ?!L9 ,	M&6M&M+M09	M#MM#)__name__
__module____qualname____doc__r   r   __annotations__r   intr    r   forward_attr__getattr__r$   r   api_boundaryrm   r   r!   r   r   r   (   s    -^ S3-S#X%% IK#s(#8@  ,,+c" S$ Z: Z:r!   r   c                     t        j                  |       D ]3  }t        |t        j                        sd}t        j                  |       y )Na  Found a JAX Tracer object passed as an argument to a custom_dce function in a position indicated by static_argnums as static. Tracers cannot be passed as static arguments to custom_dce functions; instead, static_argnums should only be used for arguments that can't be or contain JAX tracers, e.g. function-valued arguments. In particular, array-valued arguments should typically not be indicated as static_argnums.)r   tree_leaves
isinstancer   Tracerr	   UnexpectedTracerError)rh   leafmsgs      r   rM   rM      sH    ##A& .d$$A 
 ((--.r!   f.r8   r>   rC   r*   rA   c                 |   t        j                  ||      }	t        j                  ||      }
 | |	g|
 }t        |t              r5t	        |      t	        t        j
                  |            k(  rt        |      }t               t        j                  |t               g|j                  z        }t        j                  t        j                  |      d         \  }}g fd}t        j                  |      \  }}	 t        j                  |||d        g }t!        |||      D ]  \  }}}}|s|u r9t        d| d| dt        j"                  |       d|j%                          d	      t'        j(                  |t'        j*                  |      x}      sMt        d| d| d| dt        j"                  |       d|j%                          d	|j%                          d      |j-                  |        |j/                  |       |S # t        $ r  t        d| d| d| d| d	| d
| d      d w xY w)Nr   c                     t        t        j                  |      d         }| |j                  g|z         | S | j                  | g|z         | S )Nr   )r5   r   rS   extend)rh   d
num_leavesrl   sentinels      r   appendz flatten_dce_rule.<locals>.append  s_    Y++A.q12JyQ]ooxj:-. H 
ooqcJ&'Hr!   c                 
    | d u S r   r   )rh   s    r   <lambda>z"flatten_dce_rule.<locals>.<lambda>  s
    T	 r!   )is_leafzCustom DCE rule z for function za must produce an output with the same container (pytree) structure as the output of the function z'. But, the DCE rule returned structure z, while z
 returned .zy must produce values for all of the required outputs (as specified by the used_outs argument to the rule). But, at outputzC, the DCE rule returned None instead of an output with shape/dtype zR must produce an output with the same shapes/dtypes as the output of the function z. But, at outputz3, the DCE rule returned an output with shape/dtype z was expected.)r   r\   ry   listr5   treedef_childrenrX   objectr   r   unzip2tree_flatten_with_pathrS   tree_map
ValueErrorrJ   zipkeystr	str_shortr   	typematchrU   r   r8   )r~   r8   r>   rC   r*   r@   rB   rA   rf   py_used_outspy_argspy_outdummykeypathsr;   r   	rule_treeresultskpusedavalvalaval_rl   r   s                          @@r   r.   r.      sk    ))(I>,$$Wi8'\$G$&#f+  *2 # 6]F X(

"
"8fhZ(:M:M-M
N%I<<UCAFG+(A( ''/,!Yvvu6IJ ' 9iJ b$c
hYK~hZ @ ##B'( )++/>>+;*<A	?  >>$s); ;<YK~hZ @QJ&y'7'7';&< =1161B0C8^^n	.  NN3'* ++i	.? 
 

9+^H: > z!H+XhZz(1	>
 s   H )H;rG   c                 2     t        j                  |       | S r   )r   jaxpr_as_funrG   r]   r;   s      r   custom_dce_implr   2  s    	%		9	%t	,,r!   c                 4    ~| j                   | j                  fS r   )rA   effectsr   s      r   custom_dce_abstract_evalr   6  s    
			i//	//r!   	axis_datarF   rD   c          
           |D cg c]  }|t         j                  u c}t        ||      D cg c]!  \  }}}|rt        j                  ||d      n|# }}}}t        j                  | d      \  }	t
        j                  dt        dt        t        j                  t        t           f   f fd       }
t        j                  ||	|
d}D cg c]  }|rdnt         j                   }}||fS c c}w c c}}}w c c}w )Nr   Fr*   r"   c            
      .    	|  \  }}t        |       D cg c]
  \  }}|s	| }}}t        j                  |t        |
d        D cg c]
  \  }}|s	| c}}|      \  }}t        d t        ||      D              sJ ||fS c c}}w c c}}w )Nc              3   ,   K   | ]  \  }}||k(    y wr   r   ).0abs      r   	<genexpr>zGcustom_dce_batching.<locals>.batched_dce_jaxpr_thunk.<locals>.<genexpr>`  s     I$!QqAvIs   )r   r   batch_jaxprall)r*   r4   r=   r   r   used_out_batcheddce_jaxpr_batcheddce_out_batchedr   rD   
in_batchedrF   out_batcheds           r   batched_dce_jaxpr_thunkz4custom_dce_batching.<locals>.batched_dce_jaxpr_thunkM  s     *95Ix),Y)DMgdAMM)1)=)=h
:;(?@IwtQDI	*& I#o7G"HIIIIh&& N 	Js   
BB
BBrE   )r   
not_mappedr   moveaxisr   r/   rV   rW   rX   r   r7   r   r2   r[   )r   r]   dimsrF   rG   rD   r   r   rh   batched_fun_jaxprr   rl   out_dimsr   r   s   `  ` `       @@r   custom_dce_batchingr   ;  s)    7;;,,,;* T40
 

!Q %&h1a 1,
$ 
 $,#7#7J$ [ ;;''Tx~-.' ' '* !-	( 8CC!1a(---C(C	8	M <
H Ds   C?&DDc                   |D cg c]  }t        |t        j                          }}t        ||      D cg c]
  \  }}|s	| }}}t        j                  ||d      \  }}t        j                  j                  t        j                  t        j                  |      |j                  j                        g| | }	t        j                  |	t        |      g      \  }
}t!        |      }t        |
|      D cg c]1  \  }}|rt#        |      nt        j                  j%                  |      3 }}}|
|fS c c}w c c}}w c c}}w )NFr'   )ry   r   Zeror   	jvp_jaxprr   call_pr[   rP   rQ   r   ri   r(   r   
split_listr5   iternextfrom_primal_value)primalstangentsrG   r;   tin_nznzr   out_nzoutout_primalsout_tangentsout_tangents_iterps                 r   custom_dce_jvpr   m  s.   /7
8!z!RWW%%
8%
8 18EBRa8(8ll9eU;)V 	ll4$$Y/'oo88:	<C	FN	#
 #oocCK=A+|<( {F+
!R "$d)B)B1)EE,  
l	""+ 98 s   "D<
EE?6Eeqnc           	      x    t               s0t        j                  |      sdgt        |j                        z  d fS t               rdgt        |j                        z  |fS |j                  d   }|j                  d      \  t        j                  dt        dt        t        j                  t        t           f   f fd       }t        j                  |j                  |g      \  }}t        |      D cg c]
  \  }}|s	| }}}t         |j                         D cg c]
  \  }}|s	| }}}t#        |j                  d|	      }	t        j$                  ||t&        |	j(                  |j*                  |j,                        }
dg|j                  d   z  z   |
fS c c}}w c c}}w )
NFTrF   rD   new_used_outsr"   c                      t        |       rfS t        j                  	dgt        	      t        |       z
  z  |       } | \  }}t        j                  |      \  }}t        |      rJ ||fS )NF)r   r   merge_listsr5   partition_listany)
r   all_used_outs	new_jaxprall_used_insnot_usednew_used_insrD   ri   r=   r*   s
         r   new_dce_jaxpr_thunkz,custom_dce_rule.<locals>.new_dce_jaxpr_thunk  s     =H_$$	3y>C$667M
 .}=I|!00<HHl8}l""r!   r   rE   )r   r/   has_effectsr5   invarsr   paramsrV   rW   rX   r   r7   r   r   r   r   r6   dictnew_jaxpr_eqnr2   r   source_infoctx)r*   r   rF   r   r;   r   r   vr6   
new_paramsnew_eqnrD   ri   r=   s   `          @@@r   r&   r&     s   	Ys 37S_$d**^6C

O#S((zz,'*JJ01/#Y//%;;##Tx~-.# # oocjj:,7)!V 62;'$dA;&;!)S[[9B74TQB'B	jj)	* mm	oo	gg' 3::l+	+h	6	??# <Bs   
F0F00
F6;F6custom_dce_callT)multiple_resultsri   oldnewr"   c                     g }| j                   D ]C  }|j                  |u r"|j                  |j                  |             3|j                  |       E | j                  |      S )N)	primitive)eqns)r   r   r   replace)ri   r   r   new_eqnsr   s        r   r1   r1     s^     (ZZ c
}}oockkCk01ooc	
 
H	%%r!   dce_sential)?collections.abcr   r   r   typingr   jax._srcr   r   r   r	   r
   rP   r   r   r   r   jax._src.interpretersr   r   r   r   r/   register_exclusion__file__safe_mapmap
unsafe_mapsafe_zipr   
unsafe_zipregister_custom_decorator_typer   rM   transformation_with_aux2StorestrrW   AbstractValuer.   r7   r   r   AxisDatars   rX   r   r   JaxprEqnr&   	Primitiver2   r   def_impldef_effectful_abstract_evalregister_lowering	lower_funfancy_primitive_batchersprimitive_jvps	dce_rulesJaxprr1   r3   r   r!   r   <module>r     s   /     $  & % #   $ * & 4 #  # #H - ! ! !( +--Z--Z //k: k: 0k:\.  GSG88G G 	G
 ~G **+G GT-d&6&6 -0t/?/? 0
/  /
 / / c5)9)98D>)I#JJK/d#D4D4D #2,@x~ ,@DMM ,@^ t~~/0 $    o &  ( ()A B   .$..4H 3F ! !, /"0  , ,\ 	&::	& NN	&15	&	ZZ	& }-!%     '  ) )*B Cr!   