From 9a163bb4d8320f1842e9de436b719a4cc7c42bff Mon Sep 17 00:00:00 2001 From: Alessandro Fasse Date: Fri, 2 May 2025 14:46:33 +0200 Subject: [PATCH 01/24] added steps=n logic --- diffrax/_integrate.py | 82 ++++++++++++++++++------- diffrax/_saveat.py | 45 ++++++++------ test/test_event.py | 4 +- test/test_integrate.py | 1 + test/test_saveat_solution.py | 115 +++++++++++++++++++++++++++++++++-- 5 files changed, 203 insertions(+), 44 deletions(-) diff --git a/diffrax/_integrate.py b/diffrax/_integrate.py index 17461ba4..029c07ac 100644 --- a/diffrax/_integrate.py +++ b/diffrax/_integrate.py @@ -236,6 +236,7 @@ def _save( fn: Callable, save_state: SaveState, repeat: int, + pred=True, ) -> SaveState: ts = save_state.ts ys = save_state.ys @@ -244,12 +245,17 @@ def _save( ts = lax.dynamic_update_slice_in_dim( ts, jnp.broadcast_to(t, (repeat,)), save_index, axis=0 ) + y_to_save = lax.cond( + pred, + lambda: fn(t, y, args), + lambda: jtu.tree_map(lambda ys_: ys_[save_index], ys), + ) ys = jtu.tree_map( lambda ys_, y_: lax.dynamic_update_slice_in_dim( ys_, jnp.broadcast_to(y_, (repeat, *y_.shape)), save_index, axis=0 ), ys, - fn(t, y, args), + y_to_save, ) save_index = save_index + repeat @@ -482,14 +488,32 @@ def maybe_inplace(i, u, x): return eqxi.buffer_at_set(x, i, u, pred=keep_step) def save_steps(subsaveat: SubSaveAt, save_state: SaveState) -> SaveState: - if subsaveat.steps: + if subsaveat.steps != 0: + save_step = (state.num_accepted_steps % subsaveat.steps) == 0 + should_save = keep_step & save_step + + def save_fn(tprev, y, args): + return subsaveat.fn(tprev, y, args) + # TODO: Enable this, but I am not sure if possible? How do we know + # the output shape of `.fn`? We should do a dummy call to it? + if subsaveat.steps == 1: + return subsaveat.fn(tprev, y, args) + else: + return lax.cond( + should_save, + lambda: subsaveat.fn(tprev, y, args), + lambda: jtu.tree_map( + lambda y: jnp.zeros(y.shape[1:], y.dtype), save_state.ys + ), + ) + ts = maybe_inplace(save_state.save_index, tprev, save_state.ts) ys = jtu.tree_map( ft.partial(maybe_inplace, save_state.save_index), - subsaveat.fn(tprev, y, args), + save_fn(tprev, y, args), save_state.ys, ) - save_index = save_state.save_index + jnp.where(keep_step, 1, 0) + save_index = save_state.save_index + jnp.where(should_save, 1, 0) save_state = eqx.tree_at( lambda s: [s.ts, s.ys, s.save_index], save_state, @@ -500,7 +524,6 @@ def save_steps(subsaveat: SubSaveAt, save_state: SaveState) -> SaveState: save_state = jtu.tree_map( save_steps, saveat.subs, save_state, is_leaf=_is_subsaveat ) - if saveat.dense: dense_ts = maybe_inplace(dense_save_index + 1, tprev, dense_ts) dense_infos = jtu.tree_map( @@ -800,20 +823,33 @@ def _save_if_t0_equals_t1(subsaveat: SubSaveAt, save_state: SaveState) -> SaveSt ) def _save_t1(subsaveat, save_state): + if subsaveat.steps == 0: + # We're not saving the final value via `steps`, + # so we might need to save it via `t1`. + t1_saved_via_steps = False + elif subsaveat.steps == 1: + # We're definitely saving the final value via `steps`, + # so we can skip saving it via `t1`. + t1_saved_via_steps = True + else: + # We might be saving the final value via `steps`, + # so we might need to save it via `t1`. + t1_saved_via_steps = final_state.num_accepted_steps % subsaveat.steps == 0 if event is None or event.root_finder is None: - if subsaveat.t1 and not subsaveat.steps: - # If subsaveat.steps then the final value is already saved. - save_state = _save( - tfinal, yfinal, args, subsaveat.fn, save_state, repeat=1 - ) + if type(t1_saved_via_steps) is bool: + t1_not_saved_via_steps = not t1_saved_via_steps + else: + t1_not_saved_via_steps = jnp.logical_not(t1_saved_via_steps) + pred = subsaveat.t1 & t1_not_saved_via_steps else: - if subsaveat.t1 or subsaveat.steps: - # In this branch we need to replace the last value with tfinal - # and yfinal returned by the root finder also if subsaveat.steps - # because we deleted the last value after the event time above. - save_state = _save( - tfinal, yfinal, args, subsaveat.fn, save_state, repeat=1 - ) + # If we're using an event with a root finder, and are saving steps, + # then we need to write the final value here because we deleted the + # last value after the event time above. + pred = subsaveat.t1 | t1_saved_via_steps + if pred is not False: + save_state = _save( + tfinal, yfinal, args, subsaveat.fn, save_state, repeat=1, pred=pred + ) return save_state save_state = jtu.tree_map(_save_t1, saveat.subs, save_state, is_leaf=_is_subsaveat) @@ -1215,16 +1251,20 @@ def _allocate_output(subsaveat: SubSaveAt) -> SaveState: out_size += 1 if subsaveat.ts is not None: out_size += len(subsaveat.ts) - if subsaveat.steps: + if subsaveat.steps != 0: # We have no way of knowing how many steps we'll actually end up taking, and # XLA doesn't support dynamic shapes. So we just have to allocate the # maximum amount of steps we can possibly take. if max_steps is None: raise ValueError( - "`max_steps=None` is incompatible with saving at `steps=True`" + "`max_steps=None` is incompatible with saving at `steps=n`" ) - out_size += max_steps - if subsaveat.t1 and not subsaveat.steps: + out_size += max_steps // subsaveat.steps + if subsaveat.t1 and ( + (max_steps is None) + or (subsaveat.steps == 0) + or (max_steps % subsaveat.steps != 0) + ): out_size += 1 saveat_ts_index = 0 save_index = 0 diff --git a/diffrax/_saveat.py b/diffrax/_saveat.py index e0786593..d32e2b33 100644 --- a/diffrax/_saveat.py +++ b/diffrax/_saveat.py @@ -11,15 +11,6 @@ def save_y(t, y, args): return y -def _convert_ts( - ts: None | Sequence[RealScalarLike] | Real[Array, " times"], -) -> Real[Array, " times"] | None: - if ts is None or len(ts) == 0: - return None - else: - return jnp.asarray(ts) - - class SubSaveAt(eqx.Module): """Used for finer-grained control over what is saved. A PyTree of these should be passed to `SaveAt(subs=...)`. @@ -28,11 +19,29 @@ class SubSaveAt(eqx.Module): relatively niche feature and most users will probably not need to use `SubSaveAt`.) """ - t0: bool = False - t1: bool = False - ts: Real[Array, " times"] | None = eqx.field(default=None, converter=_convert_ts) - steps: bool = False - fn: Callable = save_y + t0: bool + t1: bool + ts: Real[Array, " times"] | None + steps: int + fn: Callable + + def __init__( + self, + *, + t0: bool = False, + t1: bool = False, + ts: None | Sequence[RealScalarLike] | Real[Array, " times"] = None, + steps: bool | int = 0, + fn: Callable = save_y, + ): + self.t0 = t0 + self.t1 = t1 + self.ts = jnp.asarray(ts) if ts is not None and len(ts) > 0 else None + if isinstance(steps, bool): + self.steps = 1 if steps else 0 + else: + self.steps = steps + self.fn = fn def __check_init__(self): if not self.t0 and not self.t1 and self.ts is None and not self.steps: @@ -44,7 +53,8 @@ def __check_init__(self): - `t0`: If `True`, save the initial input `y0`. - `t1`: If `True`, save the output at `t1`. - `ts`: Some array of times at which to save the output. -- `steps`: If `True`, save the output at every step of the numerical solver. +- `steps`: If `n>0`, save the output at every `n`th step of the numerical solver. + `0` means no saving. - `fn`: A function `fn(t, y, args)` which specifies what to save into `sol.ys` when using `t0`, `t1`, `ts` or `steps`. Defaults to `fn(t, y, args) -> y`, so that the evolving solution is saved. This can be useful to save only statistics of your @@ -71,7 +81,7 @@ def __init__( t0: bool = False, t1: bool = False, ts: None | Sequence[RealScalarLike] | Real[Array, " times"] = None, - steps: bool = False, + steps: bool | int = False, fn: Callable = save_y, subs: PyTree[SubSaveAt] = None, dense: bool = False, @@ -100,7 +110,8 @@ def __init__( - `t0`: If `True`, save the initial input `y0`. - `t1`: If `True`, save the output at `t1`. - `ts`: Some array of times at which to save the output. -- `steps`: If `True`, save the output at every step of the numerical solver. +- `steps`: If `n>0`, save the output at every `n`th step of the numerical solver. + `0` means no saving. - `dense`: If `True`, save dense output, that can later be evaluated at any part of the interval $[t_0, t_1]$ via `sol = diffeqsolve(...); sol.evaluate(...)`. diff --git a/test/test_event.py b/test/test_event.py index 80f0102c..12581c53 100644 --- a/test/test_event.py +++ b/test/test_event.py @@ -564,7 +564,7 @@ def cond_fn_2(t, y, args, **kwargs): @pytest.mark.parametrize("steps", (1, 2, 3, 4, 5)) -def test_event_save_steps(steps): +def test_event_save_all_steps(steps): term = diffrax.ODETerm(lambda t, y, args: (1.0, 1.0)) solver = diffrax.Tsit5() t0 = 0 @@ -604,7 +604,7 @@ def run(saveat): num_steps = [steps, steps, steps + 1, steps] yevents = [(thr, 0), (thr, 0), (thr, 0), (thr, thr)] - for saveat, n, yevent in zip(saveats, num_steps, yevents): + for saveat, n, yevent in zip(saveats, num_steps, yevents, strict=True): ts, ys = run(saveat) xs, zs = ys xevent, zevent = yevent diff --git a/test/test_integrate.py b/test/test_integrate.py index 15d83f3e..cfcaadfd 100644 --- a/test/test_integrate.py +++ b/test/test_integrate.py @@ -334,6 +334,7 @@ def get_dt_and_controller(level): diffrax.SaveAt(t1=True), diffrax.SaveAt(ts=[3.5, 0.7]), diffrax.SaveAt(steps=True), + diffrax.SaveAt(steps=2), diffrax.SaveAt(dense=True), ), ) diff --git a/test/test_saveat_solution.py b/test/test_saveat_solution.py index 8ddca38d..0f87b090 100644 --- a/test/test_saveat_solution.py +++ b/test/test_saveat_solution.py @@ -111,7 +111,7 @@ def test_saveat_solution(): assert sol.stats["num_steps"] > 0 assert sol.result == diffrax.RESULTS.successful - saveat = diffrax.SaveAt(steps=True) + saveat = diffrax.SaveAt(steps=1) sol = _integrate(saveat) assert sol.t0 == _t0 assert sol.t1 == _t1 @@ -131,6 +131,49 @@ def test_saveat_solution(): assert sol.stats["num_steps"] > 0 assert sol.result == diffrax.RESULTS.successful + saveat = diffrax.SaveAt(steps=2) + sol = _integrate(saveat) + assert sol.t0 == _t0 + assert sol.t1 == _t1 + n = (4096 - 1) // 2 + 1 + assert sol.ts.shape == (n,) # pyright: ignore + assert sol.ys.shape == (n, 1) # pyright: ignore + _ts = jnp.where(sol.ts == jnp.inf, jnp.nan, sol.ts) + with jax.numpy_rank_promotion("allow"): + _ys = _y0 * jnp.exp(-0.5 * (_ts - _t0))[:, None] + _ys = jnp.where(jnp.isnan(_ys), jnp.inf, _ys) + assert tree_allclose(sol.ys, _ys) + assert sol.controller_state is None + assert sol.solver_state is None + with pytest.raises(ValueError): + sol.evaluate(0.2, 0.8) + with pytest.raises(ValueError): + sol.derivative(0.2) + assert sol.stats["num_steps"] > 0 + assert sol.result == diffrax.RESULTS.successful + + saveat = diffrax.SaveAt(steps=2, t1=True) + sol = _integrate(saveat) + assert sol.t0 == _t0 + assert sol.t1 == _t1 + n = (4096 - 1) // 2 + 1 + assert sol.ts.shape == (n,) # pyright: ignore + assert sol.ys.shape == (n, 1) # pyright: ignore + _ts = jnp.where(sol.ts == jnp.inf, jnp.nan, sol.ts) + with jax.numpy_rank_promotion("allow"): + _ys = _y0 * jnp.exp(-0.5 * (_ts - _t0))[:, None] + _ys = jnp.where(jnp.isnan(_ys), jnp.inf, _ys) + print(_ys) + assert tree_allclose(sol.ys, _ys) + assert sol.controller_state is None + assert sol.solver_state is None + with pytest.raises(ValueError): + sol.evaluate(0.2, 0.8) + with pytest.raises(ValueError): + sol.derivative(0.2) + assert sol.stats["num_steps"] > 0 + assert sol.result == diffrax.RESULTS.successful + saveat = diffrax.SaveAt(dense=True) sol = _integrate(saveat) assert sol.t0 == _t0 @@ -147,6 +190,70 @@ def test_saveat_solution(): assert sol.result == diffrax.RESULTS.successful +def test_saveat_solution_skip_steps(): + def _step_integrate(saveat: diffrax.SaveAt): + term = diffrax.ODETerm(lambda t, y, args: -0.5 * y) + ts = jnp.array([0.0, 1.0, 2.0, 3.0, 4.0, 5.0, 6.0]) + sol_ts = diffrax.diffeqsolve( + term, + t0=ts[0], + t1=ts[-1], + y0=jnp.array([1.0]), + dt0=None, + solver=diffrax.Euler(), + saveat=saveat, + stepsize_controller=diffrax.StepTo(ts=ts), + max_steps=10, + ).ts + assert sol_ts is not None + return sol_ts[jnp.isfinite(sol_ts)] + + saveat = diffrax.SaveAt(steps=2) + ts = _step_integrate(saveat) + assert jnp.allclose(ts, jnp.array([1.0, 3.0, 5.0])) + saveat = diffrax.SaveAt(steps=2, t1=True) + ts = _step_integrate(saveat) + assert jnp.allclose(ts, jnp.array([1.0, 3.0, 5.0, 6.0])) + saveat = diffrax.SaveAt(steps=2, t1=True, t0=True) + ts = _step_integrate(saveat) + assert jnp.allclose(ts, jnp.array([0.0, 1.0, 3.0, 5.0, 6.0])) + saveat = diffrax.SaveAt(steps=3) + ts = _step_integrate(saveat) + assert jnp.allclose(ts, jnp.array([1.0, 4.0])) + saveat = diffrax.SaveAt(steps=3, t1=True) + ts = _step_integrate(saveat) + assert jnp.allclose(ts, jnp.array([1.0, 4.0, 6.0])) + saveat = diffrax.SaveAt(steps=3, t1=True, t0=True) + ts = _step_integrate(saveat) + assert jnp.allclose(ts, jnp.array([0.0, 1.0, 4.0, 6.0])) + + +def test_saveat_solution_skip_vs_saveat(): + ts = jnp.array([0.0, 1.0, 2.0, 3.0, 4.0, 5.0, 6.0]) + n = 2 + saveat_skip = diffrax.SaveAt(steps=n) + saveat = diffrax.SaveAt(ts=ts[::n]) + term = diffrax.ODETerm(lambda t, y, args: -0.5 * y) + + def solve(saveat): + return diffrax.diffeqsolve( + term, + t0=ts[0], + t1=ts[-1], + y0=jnp.array([1.0]), + dt0=None, + solver=diffrax.Euler(), + saveat=saveat, + stepsize_controller=diffrax.StepTo(ts=ts), + max_steps=10, + ) + + sol_skip = solve(saveat_skip) + sol = solve(saveat) + assert sol_skip.ts == sol.ts + assert sol_skip.ys == sol.ys + + @pytest.mark.parametrize("subs", [True, False]) def test_t0_eq_t1(subs): y0 = jnp.array([2.0]) @@ -164,7 +271,7 @@ def test_t0_eq_t1(subs): get2 = diffrax.SubSaveAt( t0=True, ts=ts, - steps=True, + steps=1, ) subs = (get0, get1, get2) saveat = diffrax.SaveAt(subs=subs) @@ -220,7 +327,7 @@ def _solve(tf): get2 = diffrax.SubSaveAt( t0=True, ts=ts, - steps=True, + steps=1, fn=lambda t, y, args: jnp.where(jnp.isinf(y), 3.0, 4.0), ) subs = (get0, get1, get2) @@ -294,7 +401,7 @@ def test_subsaveat(adjoint, multi_subs, with_fn, getkey): subsaveat_kwargs: dict = dict() get2 = diffrax.SubSaveAt(t0=True, ts=jnp.linspace(0.5, 1.5, 3), **subsaveat_kwargs) if multi_subs: - get0 = diffrax.SubSaveAt(steps=True, fn=lambda _, y, __: y[0]) + get0 = diffrax.SubSaveAt(steps=1, fn=lambda _, y, __: y[0]) get1 = diffrax.SubSaveAt( ts=jnp.linspace(0, 1, 5), t1=True, fn=lambda _, y, __: y[1] ) From 9873e5c57424ea099af2ca3bc4a429c8454244e3 Mon Sep 17 00:00:00 2001 From: Patrick Kidger <33688385+patrick-kidger@users.noreply.github.com> Date: Sat, 7 Jun 2025 14:46:15 +0200 Subject: [PATCH 02/24] Adjust save-every-step logic. --- diffrax/_integrate.py | 51 +++++++++++++++++----------------- test/test_saveat_solution.py | 54 +++++++++++++++++++++++------------- 2 files changed, 59 insertions(+), 46 deletions(-) diff --git a/diffrax/_integrate.py b/diffrax/_integrate.py index 029c07ac..e78c0004 100644 --- a/diffrax/_integrate.py +++ b/diffrax/_integrate.py @@ -242,9 +242,8 @@ def _save( ys = save_state.ys save_index = save_state.save_index - ts = lax.dynamic_update_slice_in_dim( - ts, jnp.broadcast_to(t, (repeat,)), save_index, axis=0 - ) + t_to_save = jnp.broadcast_to(static_select(pred, t, ts[save_index]), (repeat,)) + ts = lax.dynamic_update_slice_in_dim(ts, t_to_save, save_index, axis=0) y_to_save = lax.cond( pred, lambda: fn(t, y, args), @@ -484,33 +483,29 @@ def _body_fun(_save_state): save_ts, saveat.subs, save_state, is_leaf=_is_subsaveat ) - def maybe_inplace(i, u, x): - return eqxi.buffer_at_set(x, i, u, pred=keep_step) - def save_steps(subsaveat: SubSaveAt, save_state: SaveState) -> SaveState: if subsaveat.steps != 0: - save_step = (state.num_accepted_steps % subsaveat.steps) == 0 + save_step = (num_accepted_steps % subsaveat.steps) == 0 should_save = keep_step & save_step - def save_fn(tprev, y, args): - return subsaveat.fn(tprev, y, args) - # TODO: Enable this, but I am not sure if possible? How do we know - # the output shape of `.fn`? We should do a dummy call to it? - if subsaveat.steps == 1: - return subsaveat.fn(tprev, y, args) - else: - return lax.cond( - should_save, - lambda: subsaveat.fn(tprev, y, args), - lambda: jtu.tree_map( - lambda y: jnp.zeros(y.shape[1:], y.dtype), save_state.ys - ), - ) + if subsaveat.steps == 1: + y_to_save = subsaveat.fn(tprev, y, args) + else: + struct = eqx.filter_eval_shape(subsaveat.fn, tprev, y, args) + y_to_save = lax.cond( + eqxi.unvmap_any(should_save), + lambda: subsaveat.fn(tprev, y, args), + lambda: jtu.tree_map(jnp.zeros_like, struct), + ) - ts = maybe_inplace(save_state.save_index, tprev, save_state.ts) + ts = eqxi.buffer_at_set( + save_state.ts, save_state.save_index, tprev, pred=should_save + ) ys = jtu.tree_map( - ft.partial(maybe_inplace, save_state.save_index), - save_fn(tprev, y, args), + lambda _y, _ys: eqxi.buffer_at_set( + _ys, save_state.save_index, _y, pred=should_save + ), + y_to_save, save_state.ys, ) save_index = save_state.save_index + jnp.where(should_save, 1, 0) @@ -525,9 +520,13 @@ def save_fn(tprev, y, args): save_steps, saveat.subs, save_state, is_leaf=_is_subsaveat ) if saveat.dense: - dense_ts = maybe_inplace(dense_save_index + 1, tprev, dense_ts) + dense_ts = eqxi.buffer_at_set( + dense_ts, dense_save_index + 1, tprev, pred=keep_step + ) dense_infos = jtu.tree_map( - ft.partial(maybe_inplace, dense_save_index), + lambda _i, _is: eqxi.buffer_at_set( + _is, dense_save_index, _i, pred=keep_step + ), dense_info, dense_infos, ) diff --git a/test/test_saveat_solution.py b/test/test_saveat_solution.py index 0f87b090..0d8ad4c1 100644 --- a/test/test_saveat_solution.py +++ b/test/test_saveat_solution.py @@ -191,9 +191,12 @@ def test_saveat_solution(): def test_saveat_solution_skip_steps(): - def _step_integrate(saveat: diffrax.SaveAt): + def _step_integrate(saveat: diffrax.SaveAt, with_7: bool): term = diffrax.ODETerm(lambda t, y, args: -0.5 * y) - ts = jnp.array([0.0, 1.0, 2.0, 3.0, 4.0, 5.0, 6.0]) + if with_7: + ts = jnp.array([0.0, 1.0, 2.0, 3.0, 4.0, 5.0, 6.0, 7.0]) + else: + ts = jnp.array([0.0, 1.0, 2.0, 3.0, 4.0, 5.0, 6.0]) sol_ts = diffrax.diffeqsolve( term, t0=ts[0], @@ -208,24 +211,35 @@ def _step_integrate(saveat: diffrax.SaveAt): assert sol_ts is not None return sol_ts[jnp.isfinite(sol_ts)] - saveat = diffrax.SaveAt(steps=2) - ts = _step_integrate(saveat) - assert jnp.allclose(ts, jnp.array([1.0, 3.0, 5.0])) - saveat = diffrax.SaveAt(steps=2, t1=True) - ts = _step_integrate(saveat) - assert jnp.allclose(ts, jnp.array([1.0, 3.0, 5.0, 6.0])) - saveat = diffrax.SaveAt(steps=2, t1=True, t0=True) - ts = _step_integrate(saveat) - assert jnp.allclose(ts, jnp.array([0.0, 1.0, 3.0, 5.0, 6.0])) - saveat = diffrax.SaveAt(steps=3) - ts = _step_integrate(saveat) - assert jnp.allclose(ts, jnp.array([1.0, 4.0])) - saveat = diffrax.SaveAt(steps=3, t1=True) - ts = _step_integrate(saveat) - assert jnp.allclose(ts, jnp.array([1.0, 4.0, 6.0])) - saveat = diffrax.SaveAt(steps=3, t1=True, t0=True) - ts = _step_integrate(saveat) - assert jnp.allclose(ts, jnp.array([0.0, 1.0, 4.0, 6.0])) + ts = _step_integrate(diffrax.SaveAt(steps=2), with_7=True) + assert jnp.allclose(ts, jnp.array([2.0, 4.0, 6.0])) + ts = _step_integrate(diffrax.SaveAt(steps=2), with_7=False) + assert jnp.allclose(ts, jnp.array([2.0, 4.0, 6.0])) + + ts = _step_integrate(diffrax.SaveAt(steps=2, t1=True), with_7=True) + assert jnp.allclose(ts, jnp.array([2.0, 4.0, 6.0, 7.0])) + ts = _step_integrate(diffrax.SaveAt(steps=2, t1=True), with_7=False) + assert jnp.allclose(ts, jnp.array([2.0, 4.0, 6.0])) + + ts = _step_integrate(diffrax.SaveAt(steps=2, t1=True, t0=True), with_7=True) + assert jnp.allclose(ts, jnp.array([0.0, 2.0, 4.0, 6.0, 7.0])) + ts = _step_integrate(diffrax.SaveAt(steps=2, t1=True, t0=True), with_7=False) + assert jnp.allclose(ts, jnp.array([0.0, 2.0, 4.0, 6.0])) + + ts = _step_integrate(diffrax.SaveAt(steps=3), with_7=True) + assert jnp.allclose(ts, jnp.array([3.0, 6.0])) + ts = _step_integrate(diffrax.SaveAt(steps=3), with_7=False) + assert jnp.allclose(ts, jnp.array([3.0, 6.0])) + + ts = _step_integrate(diffrax.SaveAt(steps=3, t1=True), with_7=True) + assert jnp.allclose(ts, jnp.array([3.0, 6.0, 7.0])) + ts = _step_integrate(diffrax.SaveAt(steps=3, t1=True), with_7=False) + assert jnp.allclose(ts, jnp.array([3.0, 6.0])) + + ts = _step_integrate(diffrax.SaveAt(steps=3, t1=True, t0=True), with_7=True) + assert jnp.allclose(ts, jnp.array([0.0, 3.0, 6.0, 7.0])) + ts = _step_integrate(diffrax.SaveAt(steps=3, t1=True, t0=True), with_7=False) + assert jnp.allclose(ts, jnp.array([0.0, 3.0, 6.0])) def test_saveat_solution_skip_vs_saveat(): From 6739e19121da58d05c833d4a4810583200e86fdc Mon Sep 17 00:00:00 2001 From: Alessandro Fasse Date: Tue, 10 Jun 2025 10:51:39 +0200 Subject: [PATCH 03/24] fixed a comparison test between ::n and skip save at --- test/test_saveat_solution.py | 8 ++++---- 1 file changed, 4 insertions(+), 4 deletions(-) diff --git a/test/test_saveat_solution.py b/test/test_saveat_solution.py index 0d8ad4c1..ea6f4d64 100644 --- a/test/test_saveat_solution.py +++ b/test/test_saveat_solution.py @@ -245,7 +245,7 @@ def _step_integrate(saveat: diffrax.SaveAt, with_7: bool): def test_saveat_solution_skip_vs_saveat(): ts = jnp.array([0.0, 1.0, 2.0, 3.0, 4.0, 5.0, 6.0]) n = 2 - saveat_skip = diffrax.SaveAt(steps=n) + saveat_skip = diffrax.SaveAt(steps=n, t0=True) saveat = diffrax.SaveAt(ts=ts[::n]) term = diffrax.ODETerm(lambda t, y, args: -0.5 * y) @@ -259,13 +259,13 @@ def solve(saveat): solver=diffrax.Euler(), saveat=saveat, stepsize_controller=diffrax.StepTo(ts=ts), - max_steps=10, + max_steps=6, ) sol_skip = solve(saveat_skip) sol = solve(saveat) - assert sol_skip.ts == sol.ts - assert sol_skip.ys == sol.ys + assert jnp.allclose(sol_skip.ts, sol.ts) + assert jnp.allclose(sol_skip.ys, sol.ys) @pytest.mark.parametrize("subs", [True, False]) From 073071749a0f08e72aaff9a77a9bea1865d711a1 Mon Sep 17 00:00:00 2001 From: Alessandro Fasse Date: Fri, 13 Jun 2025 11:45:15 +0200 Subject: [PATCH 04/24] fixing linting --- test/test_saveat_solution.py | 6 ++++++ 1 file changed, 6 insertions(+) diff --git a/test/test_saveat_solution.py b/test/test_saveat_solution.py index ea6f4d64..3b3c06c2 100644 --- a/test/test_saveat_solution.py +++ b/test/test_saveat_solution.py @@ -264,6 +264,12 @@ def solve(saveat): sol_skip = solve(saveat_skip) sol = solve(saveat) + assert sol is not None + assert sol.ts is not None + assert sol.ys is not None + assert sol_skip is not None + assert sol_skip.ts is not None + assert sol_skip.ys is not None assert jnp.allclose(sol_skip.ts, sol.ts) assert jnp.allclose(sol_skip.ys, sol.ys) From 5f89ab4a9ecd39ea6f36bd344ea5ee9add05e574 Mon Sep 17 00:00:00 2001 From: Patrick Kidger <33688385+patrick-kidger@users.noreply.github.com> Date: Wed, 18 Jun 2025 22:24:35 +0200 Subject: [PATCH 05/24] Added saveat-steps+event test --- test/test_saveat_solution.py | 42 ++++++++++++++++++++++++++++++++++++ 1 file changed, 42 insertions(+) diff --git a/test/test_saveat_solution.py b/test/test_saveat_solution.py index 3b3c06c2..3db4e7ba 100644 --- a/test/test_saveat_solution.py +++ b/test/test_saveat_solution.py @@ -5,6 +5,7 @@ import equinox as eqx import jax import jax.numpy as jnp +import optimistix as optx import pytest from .helpers import tree_allclose @@ -274,6 +275,47 @@ def solve(saveat): assert jnp.allclose(sol_skip.ys, sol.ys) +def test_saveat_steps_with_event(): + def solve(saveat): + sol = diffrax.diffeqsolve( + diffrax.ODETerm(lambda t, y, args: -0.5 * y), + t0=0, + t1=5, + y0=1.0, + dt0=1, + solver=diffrax.Euler(), + saveat=saveat, + event=diffrax.Event( + cond_fn=lambda t, y, args, **k: t - 3.5, + root_finder=optx.Newton(rtol=1e-5, atol=1e-5), + ), + max_steps=6, + ) + assert sol.result == diffrax.RESULTS.event_occurred + assert sol.ts is not None + assert sol.ys is not None + return sol.ts, sol.ys + + ts1, ys1 = solve(diffrax.SaveAt(steps=2)) + assert jnp.allclose(ts1, jnp.array([2.0, 3.5, jnp.inf])) + # Computed using Euler + # y(1) = 0.5 + # y(2) = 0.25 + # y(3) = 0.125 + # y(4) = 0.0625 + # linearly interpolate => y(3.5) = 0.09375 + assert jnp.allclose(ys1, jnp.array([0.25, 0.09375, jnp.inf])) + ts2, ys2 = solve(diffrax.SaveAt(steps=2, t1=True)) + assert jnp.allclose(ts2, jnp.array([2.0, 3.5, jnp.inf])) + assert jnp.allclose(ys2, jnp.array([0.25, 0.09375, jnp.inf])) + ts3, ys3 = solve(diffrax.SaveAt(steps=3)) + assert jnp.allclose(ts3, jnp.array([3.0, jnp.inf])) + assert jnp.allclose(ys3, jnp.array([0.125, jnp.inf])) + ts4, ys4 = solve(diffrax.SaveAt(steps=3, t1=True)) + assert jnp.allclose(ts4, jnp.array([3.0, 3.5])) + assert jnp.allclose(ys4, jnp.array([0.125, 0.09375])) + + @pytest.mark.parametrize("subs", [True, False]) def test_t0_eq_t1(subs): y0 = jnp.array([2.0]) From 012994121ec95933fd4a0388205f336315ccc705 Mon Sep 17 00:00:00 2001 From: LuggiStruggi Date: Wed, 18 Jun 2025 22:42:13 +0200 Subject: [PATCH 06/24] Introduction of bidirectional vs. unidirectional triggering of events --- diffrax/_event.py | 36 ++++++++++++++++++++- diffrax/_integrate.py | 20 +++++++++--- test/test_event.py | 74 +++++++++++++++++++++++++++++++++++++++++++ 3 files changed, 125 insertions(+), 5 deletions(-) diff --git a/diffrax/_event.py b/diffrax/_event.py index 12d11114..eff74e24 100644 --- a/diffrax/_event.py +++ b/diffrax/_event.py @@ -3,6 +3,7 @@ import equinox as eqx import optimistix as optx +from jax.tree import flatten, unflatten from jaxtyping import Array, PyTree from ._custom_types import BoolScalarLike, FloatScalarLike, RealScalarLike @@ -20,7 +21,36 @@ class Event(eqx.Module): """ cond_fn: PyTree[Callable[..., BoolScalarLike | RealScalarLike]] - root_finder: optx.AbstractRootFinder | None = None + trig_dir: PyTree[None | bool] + root_finder: optx.AbstractRootFinder | None + + def __init__( + self, + cond_fn, + root_finder: optx.AbstractRootFinder | None = None, + trig_dir: None | bool | PyTree[None | bool] = None, + ): + vals_cond, treedef_cond = flatten(cond_fn) + + if isinstance(trig_dir, bool) or trig_dir is None: + vals_trig = [trig_dir] * len(vals_cond) + treedef_trig = treedef_cond + else: + vals_trig, treedef_trig = flatten(trig_dir, is_leaf=lambda x: x is None) + + if treedef_cond != treedef_trig: + raise ValueError("Missmatch in the structure of cond_fn and trigger_dir") + + if not all(x is None or isinstance(x, bool) for x in vals_trig): + raise ValueError( + "`trig_dir` must be a None, bool or a PyTree of None | bools" + + " with the same structure as cond_fn" + ) + + trig_tree = unflatten(treedef_cond, vals_trig) + self.cond_fn = cond_fn + self.root_finder = root_finder + self.trig_dir = trig_tree Event.__init__.__doc__ = """**Arguments:** @@ -39,6 +69,10 @@ class Event(eqx.Module): [`optimistix.Newton`](https://docs.kidger.site/optimistix/api/root_find/#optimistix.Newton) would be a typical choice here. +- `trig_dir`: None or bool or PyTree of None or bool of the same shape as cond_fn, + that decides for each cond_fn if it triggers an event from a zero-cossing in both + directions (None), from an upcrossing (True) or from a downcrossing (False). + !!! Example Consider a bouncing ball dropped from some intial height $x_0$. We can model diff --git a/diffrax/_integrate.py b/diffrax/_integrate.py index e78c0004..545e6dd7 100644 --- a/diffrax/_integrate.py +++ b/diffrax/_integrate.py @@ -543,7 +543,7 @@ def save_steps(subsaveat: SubSaveAt, save_state: SaveState) -> SaveState: event_tnext = state.tnext event_dense_info = dense_info - def _outer_cond_fn(cond_fn_i, old_event_value_i): + def _outer_cond_fn(cond_fn_i, old_event_value_i, trig_dir_i): new_event_value_i = cond_fn_i( tprev, y, @@ -577,9 +577,19 @@ def _outer_cond_fn(cond_fn_i, old_event_value_i): f"{new_dtype}." ) if jnp.issubdtype(new_dtype, jnp.floating): - event_mask_i = jnp.sign(old_event_value_i) != jnp.sign( - new_event_value_i - ) + if trig_dir_i is None: + event_mask_i = jnp.sign(old_event_value_i) != jnp.sign( + new_event_value_i + ) + elif trig_dir_i: + event_mask_i = (jnp.sign(old_event_value_i) <= 0) & ( + jnp.sign(new_event_value_i) > 0 + ) + else: + event_mask_i = (jnp.sign(old_event_value_i) > 0) & ( + jnp.sign(new_event_value_i) <= 0 + ) + elif jnp.issubdtype(new_dtype, jnp.bool_): event_mask_i = new_event_value_i else: @@ -593,8 +603,10 @@ def _outer_cond_fn(cond_fn_i, old_event_value_i): _outer_cond_fn, event.cond_fn, state.event_values, + event.trig_dir, is_leaf=callable, ) + event_structure = jtu.tree_structure(event.cond_fn, is_leaf=callable) event_values, event_mask = jtu.tree_transpose( event_structure, diff --git a/test/test_event.py b/test/test_event.py index 12581c53..d4e37850 100644 --- a/test/test_event.py +++ b/test/test_event.py @@ -709,3 +709,77 @@ def save_fn(t, y, args): assert jnp.sum(jnp.isfinite(ts_2)) == steps assert jnp.all(jnp.isclose(ys_2[steps - 1], jnp.array([thr, 0]), atol=1e-5)) assert jnp.all(jnp.isclose(ys_1.y[ts_event - 1], last_save, atol=1e-5)) + + +def test_event_trig_dir(): + term = diffrax.ODETerm(lambda t, y, args: jnp.array([1.0, 1.0])) + solver = diffrax.Tsit5() + t0 = 0.0 + t1 = 10.0 + dt0 = 1.0 + y0 = jnp.array([0, 1]) + + def cond_fn0(t, y, args, **kwargs): + return y[0] - 5.0 + + def cond_fn1(t, y, args, **kwargs): + return y[1] - 5.0 + + root_finder = optx.Newton(1e-5, 1e-5, optx.rms_norm) + event = diffrax.Event((cond_fn0, cond_fn1), root_finder, (True, False)) + sol = diffrax.diffeqsolve(term, solver, t0, t1, dt0, y0, event=event) + + assert jnp.isclose(cast(Array, sol.ts)[-1], 5.0) + assert jnp.all(jnp.isclose(cast(Array, sol.ys)[-1], jnp.array([5.0, 6.0]))) + + +def test_event_trig_dir_single_true(): + term = diffrax.ODETerm(lambda t, y, args: jnp.array([1.0, 1.0])) + solver = diffrax.Tsit5() + t0 = 0.0 + t1 = 10.0 + dt0 = 1.0 + y0 = jnp.array([0, 1]) + + def cond_fn0(t, y, args, **kwargs): + return y[0] - 5.0 + + def cond_fn1(t, y, args, **kwargs): + return -(y[1] - 5.0) + + root_finder = optx.Newton(1e-5, 1e-5, optx.rms_norm) + event = diffrax.Event((cond_fn0, cond_fn1), root_finder, True) + sol = diffrax.diffeqsolve(term, solver, t0, t1, dt0, y0, event=event) + + assert jnp.isclose(cast(Array, sol.ts)[-1], 5.0) + assert jnp.all(jnp.isclose(cast(Array, sol.ys)[-1], jnp.array([5.0, 6.0]))) + + +def test_event_trig_dir_single_none(): + term = diffrax.ODETerm(lambda t, y, args: jnp.array([1.0, 1.0])) + solver = diffrax.Tsit5() + t0 = 0.0 + t1 = 10.0 + dt0 = 1.0 + y0 = jnp.array([0, 1]) + + def cond_fn0(t, y, args, **kwargs): + return y[0] - 5.0 + + def cond_fn1(t, y, args, **kwargs): + return -(y[1] - 5.0) + + root_finder = optx.Newton(1e-5, 1e-5, optx.rms_norm) + event = diffrax.Event((cond_fn0, cond_fn1), root_finder, None) + sol = diffrax.diffeqsolve(term, solver, t0, t1, dt0, y0, event=event) + + assert jnp.isclose(cast(Array, sol.ts)[-1], 4.0) + assert jnp.all(jnp.isclose(cast(Array, sol.ys)[-1], jnp.array([4.0, 5.0]))) + + +def test_event_trig_dir_pytree_structure(): + f = lambda x: x + root_finder = optx.Newton(1e-5, 1e-5, optx.rms_norm) + diffrax.Event([f, f, [f, (f)]], root_finder, [True, None, [False, (True)]]) + with pytest.raises(ValueError): + diffrax.Event([f, f, [f, f, f]], root_finder, [True, None, [False, (True)]]) From 1b3daf0fe1cf2b6c25f773a28672ca4ef658e5b0 Mon Sep 17 00:00:00 2001 From: Patrick Kidger <33688385+patrick-kidger@users.noreply.github.com> Date: Wed, 18 Jun 2025 22:52:34 +0200 Subject: [PATCH 07/24] Extend event crossing tests --- test/test_event.py | 101 +++++++++++++++++++++------------------------ 1 file changed, 47 insertions(+), 54 deletions(-) diff --git a/test/test_event.py b/test/test_event.py index d4e37850..ea368141 100644 --- a/test/test_event.py +++ b/test/test_event.py @@ -719,67 +719,60 @@ def test_event_trig_dir(): dt0 = 1.0 y0 = jnp.array([0, 1]) - def cond_fn0(t, y, args, **kwargs): - return y[0] - 5.0 - - def cond_fn1(t, y, args, **kwargs): - return y[1] - 5.0 - - root_finder = optx.Newton(1e-5, 1e-5, optx.rms_norm) - event = diffrax.Event((cond_fn0, cond_fn1), root_finder, (True, False)) - sol = diffrax.diffeqsolve(term, solver, t0, t1, dt0, y0, event=event) - - assert jnp.isclose(cast(Array, sol.ts)[-1], 5.0) - assert jnp.all(jnp.isclose(cast(Array, sol.ys)[-1], jnp.array([5.0, 6.0]))) - - -def test_event_trig_dir_single_true(): - term = diffrax.ODETerm(lambda t, y, args: jnp.array([1.0, 1.0])) - solver = diffrax.Tsit5() - t0 = 0.0 - t1 = 10.0 - dt0 = 1.0 - y0 = jnp.array([0, 1]) - - def cond_fn0(t, y, args, **kwargs): - return y[0] - 5.0 - - def cond_fn1(t, y, args, **kwargs): - return -(y[1] - 5.0) - - root_finder = optx.Newton(1e-5, 1e-5, optx.rms_norm) - event = diffrax.Event((cond_fn0, cond_fn1), root_finder, True) - sol = diffrax.diffeqsolve(term, solver, t0, t1, dt0, y0, event=event) - - assert jnp.isclose(cast(Array, sol.ts)[-1], 5.0) - assert jnp.all(jnp.isclose(cast(Array, sol.ys)[-1], jnp.array([5.0, 6.0]))) - - -def test_event_trig_dir_single_none(): - term = diffrax.ODETerm(lambda t, y, args: jnp.array([1.0, 1.0])) - solver = diffrax.Tsit5() - t0 = 0.0 - t1 = 10.0 - dt0 = 1.0 - y0 = jnp.array([0, 1]) - - def cond_fn0(t, y, args, **kwargs): - return y[0] - 5.0 + def up_cond(t, y, args, **kwargs): + del t, args, kwargs + y0, _ = y + return y0 - 5.0 - def cond_fn1(t, y, args, **kwargs): - return -(y[1] - 5.0) + def down_cond(t, y, args, **kwargs): + del t, args, kwargs + _, y1 = y + return -(y1 - 5.0) root_finder = optx.Newton(1e-5, 1e-5, optx.rms_norm) - event = diffrax.Event((cond_fn0, cond_fn1), root_finder, None) - sol = diffrax.diffeqsolve(term, solver, t0, t1, dt0, y0, event=event) - assert jnp.isclose(cast(Array, sol.ts)[-1], 4.0) - assert jnp.all(jnp.isclose(cast(Array, sol.ys)[-1], jnp.array([4.0, 5.0]))) + def run(event): + sol = diffrax.diffeqsolve(term, solver, t0, t1, dt0, y0, event=event) + assert sol.ts is not None + assert sol.ys is not None + [t_final] = sol.ts + [y_final] = sol.ys + return t_final, y_final + + event = diffrax.Event((up_cond, down_cond), root_finder, True) + t_final, y_final = run(event) + assert jnp.allclose(t_final, 5.0) + assert jnp.allclose(y_final, jnp.array([5.0, 6.0])) + + event = diffrax.Event((up_cond, down_cond), root_finder, False) + t_final, y_final = run(event) + assert jnp.allclose(t_final, 4.0) + assert jnp.allclose(y_final, jnp.array([4.0, 5.0])) + + event = diffrax.Event((up_cond, down_cond), root_finder, (True, True)) + t_final, y_final = run(event) + assert jnp.allclose(t_final, 5.0) + assert jnp.allclose(y_final, jnp.array([5.0, 6.0])) + + event = diffrax.Event((up_cond, down_cond), root_finder, (True, False)) + t_final, y_final = run(event) + assert jnp.allclose(t_final, 4.0) + assert jnp.allclose(y_final, jnp.array([4.0, 5.0])) + + event = diffrax.Event((up_cond, down_cond), root_finder, (False, True)) + t_final, y_final = run(event) + assert jnp.allclose(t_final, 10.0) + assert jnp.allclose(y_final, jnp.array([10.0, 11.0])) + + event = diffrax.Event((up_cond, down_cond), root_finder, (False, None)) + t_final, y_final = run(event) + assert jnp.allclose(t_final, 4.0) + assert jnp.allclose(y_final, jnp.array([4.0, 5.0])) def test_event_trig_dir_pytree_structure(): f = lambda x: x root_finder = optx.Newton(1e-5, 1e-5, optx.rms_norm) - diffrax.Event([f, f, [f, (f)]], root_finder, [True, None, [False, (True)]]) + diffrax.Event([f, f, [f, f]], root_finder, [True, None, [False, True]]) with pytest.raises(ValueError): - diffrax.Event([f, f, [f, f, f]], root_finder, [True, None, [False, (True)]]) + diffrax.Event([f, f, [f, f, f]], root_finder, [True, None, [False, True]]) From 49752b84cc43ac46a642f1aac037ece2d95179e6 Mon Sep 17 00:00:00 2001 From: Patrick Kidger <33688385+patrick-kidger@users.noreply.github.com> Date: Wed, 18 Jun 2025 23:05:24 +0200 Subject: [PATCH 08/24] Fixes for pytree-valued condition functions. --- diffrax/_event.py | 39 +++++++++++++++-------------- diffrax/_integrate.py | 8 +++--- test/test_event.py | 57 +++++++++++++++++++++++++++++++++++++++++++ 3 files changed, 80 insertions(+), 24 deletions(-) diff --git a/diffrax/_event.py b/diffrax/_event.py index eff74e24..f1d1ea7e 100644 --- a/diffrax/_event.py +++ b/diffrax/_event.py @@ -2,8 +2,8 @@ from collections.abc import Callable import equinox as eqx +import jax.tree_util as jtu import optimistix as optx -from jax.tree import flatten, unflatten from jaxtyping import Array, PyTree from ._custom_types import BoolScalarLike, FloatScalarLike, RealScalarLike @@ -21,36 +21,33 @@ class Event(eqx.Module): """ cond_fn: PyTree[Callable[..., BoolScalarLike | RealScalarLike]] - trig_dir: PyTree[None | bool] + direction: PyTree[None | bool] root_finder: optx.AbstractRootFinder | None def __init__( self, cond_fn, root_finder: optx.AbstractRootFinder | None = None, - trig_dir: None | bool | PyTree[None | bool] = None, + direction: None | bool | PyTree[None | bool] = None, ): - vals_cond, treedef_cond = flatten(cond_fn) + if direction in (None, False, True): + direction = jtu.tree_map(lambda _: direction, cond_fn, is_leaf=callable) - if isinstance(trig_dir, bool) or trig_dir is None: - vals_trig = [trig_dir] * len(vals_cond) - treedef_trig = treedef_cond - else: - vals_trig, treedef_trig = flatten(trig_dir, is_leaf=lambda x: x is None) - - if treedef_cond != treedef_trig: - raise ValueError("Missmatch in the structure of cond_fn and trigger_dir") + direction_leaves, direction_structure = jtu.tree_flatten( + direction, is_leaf=lambda x: x is None + ) + if direction_structure != jtu.tree_structure(cond_fn, is_leaf=callable): + raise ValueError("Missmatch in the structure of `cond_fn` and `direction`.") - if not all(x is None or isinstance(x, bool) for x in vals_trig): + if any(x not in (None, False, True) for x in direction_leaves): raise ValueError( - "`trig_dir` must be a None, bool or a PyTree of None | bools" - + " with the same structure as cond_fn" + "`trig_dir` must be a `None`, `bool`, or a PyTree of `None | bool`s " + "with the same structure as `cond_fn`." ) - trig_tree = unflatten(treedef_cond, vals_trig) self.cond_fn = cond_fn self.root_finder = root_finder - self.trig_dir = trig_tree + self.direction = direction Event.__init__.__doc__ = """**Arguments:** @@ -69,9 +66,11 @@ def __init__( [`optimistix.Newton`](https://docs.kidger.site/optimistix/api/root_find/#optimistix.Newton) would be a typical choice here. -- `trig_dir`: None or bool or PyTree of None or bool of the same shape as cond_fn, - that decides for each cond_fn if it triggers an event from a zero-cossing in both - directions (None), from an upcrossing (True) or from a downcrossing (False). +- `direction`: `None` or `bool` or PyTree of `None | bool` of the same shape as + `cond_fn`, that decides for each `cond_fn` if it triggers an event from a + zero-cossing in both directions (`None`), from an upcrossing (`True`) or from a + downcrossing (`False`). Only needed for those `cond_fn` which return floating point + numbers; ignored for those `cond_fn` which return booleans. !!! Example diff --git a/diffrax/_integrate.py b/diffrax/_integrate.py index 545e6dd7..8d1c66b4 100644 --- a/diffrax/_integrate.py +++ b/diffrax/_integrate.py @@ -543,7 +543,7 @@ def save_steps(subsaveat: SubSaveAt, save_state: SaveState) -> SaveState: event_tnext = state.tnext event_dense_info = dense_info - def _outer_cond_fn(cond_fn_i, old_event_value_i, trig_dir_i): + def _outer_cond_fn(cond_fn_i, old_event_value_i, direction_i): new_event_value_i = cond_fn_i( tprev, y, @@ -577,11 +577,11 @@ def _outer_cond_fn(cond_fn_i, old_event_value_i, trig_dir_i): f"{new_dtype}." ) if jnp.issubdtype(new_dtype, jnp.floating): - if trig_dir_i is None: + if direction_i is None: event_mask_i = jnp.sign(old_event_value_i) != jnp.sign( new_event_value_i ) - elif trig_dir_i: + elif direction_i: event_mask_i = (jnp.sign(old_event_value_i) <= 0) & ( jnp.sign(new_event_value_i) > 0 ) @@ -603,7 +603,7 @@ def _outer_cond_fn(cond_fn_i, old_event_value_i, trig_dir_i): _outer_cond_fn, event.cond_fn, state.event_values, - event.trig_dir, + event.direction, is_leaf=callable, ) diff --git a/test/test_event.py b/test/test_event.py index ea368141..5de0ded6 100644 --- a/test/test_event.py +++ b/test/test_event.py @@ -776,3 +776,60 @@ def test_event_trig_dir_pytree_structure(): diffrax.Event([f, f, [f, f]], root_finder, [True, None, [False, True]]) with pytest.raises(ValueError): diffrax.Event([f, f, [f, f, f]], root_finder, [True, None, [False, True]]) + + +def test_event_with_pytree_valued_condition_function(): + term = diffrax.ODETerm(lambda t, y, args: jnp.array([1.0, 1.0])) + solver = diffrax.Tsit5() + t0 = 0.0 + t1 = 10.0 + dt0 = 1.0 + y0 = jnp.array([0, 1]) + + class CondFn(eqx.Module): + crossing: tuple[tuple[float]] + downcrossing: bool + + def __call__(self, t, y, args, **kwargs): + del t, args, kwargs + y0, _ = y + [[crossing]] = self.crossing + out = y0 - crossing + if self.downcrossing: + out = -out + return out + + def another_cond_fn(t, y, args, **kwargs): + del t, y, args, kwargs + return 5.0 + + root_finder = optx.Newton(1e-5, 1e-5, optx.rms_norm) + + def run(event): + sol = diffrax.diffeqsolve(term, solver, t0, t1, dt0, y0, event=event) + assert sol.ts is not None + assert sol.ys is not None + [t_final] = sol.ts + [y_final] = sol.ys + return t_final, y_final + + event = diffrax.Event( + (CondFn(((3.0,),), False), another_cond_fn), root_finder, (True, None) + ) + t_final, y_final = run(event) + assert jnp.allclose(t_final, 3.0) + assert jnp.allclose(y_final, jnp.array([3.0, 4.0])) + + event = diffrax.Event( + (CondFn(((3.0,),), False), another_cond_fn), root_finder, (None, False) + ) + t_final, y_final = run(event) + assert jnp.allclose(t_final, 3.0) + assert jnp.allclose(y_final, jnp.array([3.0, 4.0])) + + event = diffrax.Event( + (CondFn(((3.0,),), False), another_cond_fn), root_finder, (False, False) + ) + t_final, y_final = run(event) + assert jnp.allclose(t_final, 10.0) + assert jnp.allclose(y_final, jnp.array([10.0, 11.0])) From 97b656e2386152c8cdacaa30b47287cb17fd8135 Mon Sep 17 00:00:00 2001 From: Jonathan Brodrick Date: Wed, 16 Jul 2025 02:16:45 +0100 Subject: [PATCH 09/24] Improve ConstantStepSize incrementation --- diffrax/_step_size_controller/constant.py | 63 ++++++++++++++++++++--- test/test_progress_meter.py | 10 ++-- 2 files changed, 60 insertions(+), 13 deletions(-) diff --git a/diffrax/_step_size_controller/constant.py b/diffrax/_step_size_controller/constant.py index e4100f83..a7e5c879 100644 --- a/diffrax/_step_size_controller/constant.py +++ b/diffrax/_step_size_controller/constant.py @@ -11,7 +11,15 @@ from .base import AbstractStepSizeController -class ConstantStepSize(AbstractStepSizeController[RealScalarLike, RealScalarLike]): +# ConstantStepSizeState = (steps_completed, num_steps, t0_sim, t1_sim_or_dt0) +_ConstantStepSizeState = tuple[ + IntScalarLike, IntScalarLike, RealScalarLike, RealScalarLike +] + + +class ConstantStepSize( + AbstractStepSizeController[_ConstantStepSizeState, RealScalarLike] +): """Use a constant step size, equal to the `dt0` argument of [`diffrax.diffeqsolve`][]. """ @@ -29,14 +37,23 @@ def init( args: Args, func: Callable[[PyTree[AbstractTerm], RealScalarLike, Y, Args], VF], error_order: RealScalarLike | None, - ) -> tuple[RealScalarLike, RealScalarLike]: - del terms, t1, y0, args, func, error_order + ) -> tuple[RealScalarLike, _ConstantStepSizeState]: + del terms, y0, args, func, error_order if dt0 is None: raise ValueError( "Constant step size solvers cannot select step size automatically; " "please pass a value for `dt0`." ) - return t0 + dt0, dt0 + steps_completed = jnp.asarray(1, dtype=jnp.int32) + # Special case for infinite t1, allow termination based on conditional tests + # Use num_steps=-1 to ensure finite int + num_steps = jnp.where( + jnp.isfinite(t1), + jnp.astype(jnp.ceil((t1 - t0) / eqxi.nextafter(dt0)), jnp.int32), + -1, + ) + t1_sim_or_dt0 = jnp.where(jnp.isfinite(t1), t1, dt0) + return t0 + dt0, (steps_completed, num_steps, t0, t1_sim_or_dt0) def adapt_step_size( self, @@ -47,15 +64,45 @@ def adapt_step_size( args: Args, y_error: Y | None, error_order: RealScalarLike | None, - controller_state: RealScalarLike, - ) -> tuple[bool, RealScalarLike, RealScalarLike, bool, RealScalarLike, RESULTS]: + controller_state: _ConstantStepSizeState, + ) -> tuple[ + bool, + RealScalarLike, + RealScalarLike, + bool, + _ConstantStepSizeState, + RESULTS, + ]: del t0, y0, y1_candidate, args, y_error, error_order + steps_already_completed, num_steps, t0_sim, t1_sim_or_dt0 = controller_state + # Number of steps that will be completed when this function returns. + steps_completed = steps_already_completed + 1 + + time_dtype = jnp.result_type(t0_sim, t1_sim_or_dt0) + + # Calculate step size by calculating fraction of `t1 - t0` to avoid compounding + # of truncation/rounding errors + t1_next = jnp.where( + num_steps >= 0, + jnp.where( + steps_completed == num_steps, + t1_sim_or_dt0, + t0_sim + + (t1_sim_or_dt0 - t0_sim) + * jnp.astype(steps_completed, time_dtype) + / jnp.astype(num_steps, time_dtype), + ), + # Special case for non-finite t1_sim + # in this t1_sim_or_dt0 is dt0 + t1 + t1_sim_or_dt0, + ) + return ( True, t1, - t1 + controller_state, + t1_next, False, - controller_state, + (steps_completed, num_steps, t0_sim, t1_sim_or_dt0), RESULTS.successful, ) diff --git a/test/test_progress_meter.py b/test/test_progress_meter.py index a9613c9e..07db1277 100644 --- a/test/test_progress_meter.py +++ b/test/test_progress_meter.py @@ -53,31 +53,31 @@ def solve(t0): t1=5, dt0=0.01, y0=1.0, - progress_meter=diffrax.TextProgressMeter(minimum_increase=0.1), + progress_meter=diffrax.TextProgressMeter(minimum_increase=0.0999), ) solve(2.0) jax.effects_barrier() captured = capfd.readouterr() - expected = "0.00%\n10.33%\n20.67%\n31.00%\n41.33%\n51.67%\n62.00%\n72.33%\n82.67%\n93.00%\n100.00%\n" # noqa: E501 + expected = "%\n".join(f"{x:.2f}" for x in jnp.linspace(0, 100, num=11)) + "%\n" assert captured.out == expected jax.vmap(solve)(jnp.arange(3.0)) jax.effects_barrier() captured = capfd.readouterr() - expected = "0.00%\n10.00%\n20.00%\n30.00%\n40.00%\n50.20%\n60.40%\n70.60%\n80.80%\n91.00%\n100.00%\n" # noqa: E501 + expected = "%\n".join(f"{x:.2f}" for x in jnp.linspace(0, 100, num=11)) + "%\n" assert captured.out == expected jax.jit(solve)(2.0) jax.effects_barrier() captured = capfd.readouterr() - expected = "0.00%\n10.33%\n20.67%\n31.00%\n41.33%\n51.67%\n62.00%\n72.33%\n82.67%\n93.00%\n100.00%\n" # noqa: E501 + expected = "%\n".join(f"{x:.2f}" for x in jnp.linspace(0, 100, num=11)) + "%\n" assert captured.out == expected jax.jit(jax.vmap(solve))(jnp.arange(3.0)) jax.effects_barrier() captured = capfd.readouterr() - expected = "0.00%\n10.00%\n20.00%\n30.00%\n40.00%\n50.20%\n60.40%\n70.60%\n80.80%\n91.00%\n100.00%\n" # noqa: E501 + expected = "%\n".join(f"{x:.2f}" for x in jnp.linspace(0, 100, num=11)) + "%\n" assert captured.out == expected From a1305ea4f689b4e8db187c118a0eb1909ba9966b Mon Sep 17 00:00:00 2001 From: Patrick Kidger <33688385+patrick-kidger@users.noreply.github.com> Date: Wed, 30 Jul 2025 20:32:23 +0200 Subject: [PATCH 10/24] Tweaked layout of ConstantStepSize code --- diffrax/_step_size_controller/constant.py | 43 +++++++++++------------ 1 file changed, 20 insertions(+), 23 deletions(-) diff --git a/diffrax/_step_size_controller/constant.py b/diffrax/_step_size_controller/constant.py index a7e5c879..02252869 100644 --- a/diffrax/_step_size_controller/constant.py +++ b/diffrax/_step_size_controller/constant.py @@ -45,13 +45,12 @@ def init( "please pass a value for `dt0`." ) steps_completed = jnp.asarray(1, dtype=jnp.int32) - # Special case for infinite t1, allow termination based on conditional tests - # Use num_steps=-1 to ensure finite int - num_steps = jnp.where( - jnp.isfinite(t1), - jnp.astype(jnp.ceil((t1 - t0) / eqxi.nextafter(dt0)), jnp.int32), - -1, - ) + # `eqxi.nextafter` to handle floating point error, see + # https://github.com/patrick-kidger/diffrax/pull/666#discussion_r2215868590 + num_steps = jnp.astype(jnp.ceil((t1 - t0) / eqxi.nextafter(dt0)), jnp.int32) + # Use `num_steps=-1` as a marker to indicate that `diffeqsolve(..., t1=...)` is + # infinite. + num_steps = jnp.where(jnp.isfinite(t1), num_steps, -1) t1_sim_or_dt0 = jnp.where(jnp.isfinite(t1), t1, dt0) return t0 + dt0, (steps_completed, num_steps, t0, t1_sim_or_dt0) @@ -78,25 +77,23 @@ def adapt_step_size( # Number of steps that will be completed when this function returns. steps_completed = steps_already_completed + 1 + # Calculate step size by calculating fraction of `t1 - t0` -- rather than just + # adding up `dt0` multiple times -- to avoid compounding of truncation/rounding + # errors. time_dtype = jnp.result_type(t0_sim, t1_sim_or_dt0) - - # Calculate step size by calculating fraction of `t1 - t0` to avoid compounding - # of truncation/rounding errors - t1_next = jnp.where( - num_steps >= 0, - jnp.where( - steps_completed == num_steps, - t1_sim_or_dt0, - t0_sim - + (t1_sim_or_dt0 - t0_sim) - * jnp.astype(steps_completed, time_dtype) - / jnp.astype(num_steps, time_dtype), - ), - # Special case for non-finite t1_sim - # in this t1_sim_or_dt0 is dt0 - t1 + t1_sim_or_dt0, + t1_next = t0_sim + (t1_sim_or_dt0 - t0_sim) * ( + jnp.astype(steps_completed, time_dtype) / jnp.astype(num_steps, time_dtype) ) + # If we're on the final step then use `t1` directly, this time to avoid + # floating-point weirdness in the above. (Not sure if necessary?) + t1_next = jnp.where(steps_completed == num_steps, t1_sim_or_dt0, t1_next) + + # If `num_steps == -1` then we use that as a marker to indicate that we have an + # infinite `diffeqsolve(..., t1=...)`. In this case then never mind everything + # above, we really do just want to keep adding on `dt0` multiple times. + t1_next = jnp.where(num_steps >= 0, t1_next, t1 + t1_sim_or_dt0) + return ( True, t1, From 5d9f6b99af175ac3c0631b81a5af4d3de71b49ad Mon Sep 17 00:00:00 2001 From: Jonathan Brodrick Date: Sat, 2 Aug 2025 20:01:04 +0100 Subject: [PATCH 11/24] Use 100 ULP's to clip timesteps close to t1 (#660) * Use 100 ULP's to clip timesteps close to t1 * test that t1-t0 > 100 ULP's * revert testing as t1 is traced * remove unnecessary pyright ignores --- diffrax/_integrate.py | 19 ++++++++++--------- 1 file changed, 10 insertions(+), 9 deletions(-) diff --git a/diffrax/_integrate.py b/diffrax/_integrate.py index 8d1c66b4..bc319d40 100644 --- a/diffrax/_integrate.py +++ b/diffrax/_integrate.py @@ -263,15 +263,11 @@ def _save( ) -def _clip_to_end(tprev, tnext, t1, keep_step): - # The tolerance means that we don't end up with too-small intervals for - # dense output, which then gives numerically unstable answers due to floating +def _clip_to_end(tprev, tnext, t1, t1_clip_floor, keep_step): + # The tolerance of ~100 ULP's means that we don't end up with too-small intervals + # for dense output, which then gives numerically unstable answers due to floating # point errors. - if tnext.dtype == jnp.dtype("float64"): - tol = 1e-10 - else: - tol = 1e-6 - clip = tnext > t1 - tol + clip = tnext > t1_clip_floor tclip = jnp.where(keep_step, t1, tprev + 0.5 * (t1 - tprev)) return jnp.where(clip, tclip, tnext) @@ -308,6 +304,11 @@ def loop( outer_while_loop, progress_meter, ): + # Calculate in advance t1 - 100 ULP's: the threshold at which to round tnext to t1 + t1_clip_floor = t1 + for _ in range(100): + t1_clip_floor = eqxi.prevbefore(t1_clip_floor) + if saveat.dense: dense_ts = init_state.dense_ts dense_ts = dense_ts.at[0].set(t0) @@ -397,7 +398,7 @@ def body_fun_aux(state): # tprev = jnp.minimum(tprev, t1) - tnext = _clip_to_end(tprev, tnext, t1, keep_step) + tnext = _clip_to_end(tprev, tnext, t1, t1_clip_floor, keep_step) progress_meter_state = progress_meter.step( state.progress_meter_state, linear_rescale(t0, tprev, t1) From 1ac005dde8862087a05c83b82cdcb1ad608b34da Mon Sep 17 00:00:00 2001 From: Patrick Kidger <33688385+patrick-kidger@users.noreply.github.com> Date: Sun, 13 Jul 2025 12:07:23 +0200 Subject: [PATCH 12/24] Tests that a jump at t1 is saved. --- test/test_adaptive_stepsize_controller.py | 25 +++++++++++++++++++++++ 1 file changed, 25 insertions(+) diff --git a/test/test_adaptive_stepsize_controller.py b/test/test_adaptive_stepsize_controller.py index 8161b9d8..21c24e4a 100644 --- a/test/test_adaptive_stepsize_controller.py +++ b/test/test_adaptive_stepsize_controller.py @@ -336,3 +336,28 @@ def test_implicit_solver_with_clip_controller(new: bool): max_steps=16384, saveat=diffrax.SaveAt(t1=True), ) + + +# https://github.com/patrick-kidger/diffrax/issues/663 +# `jump_ts` sets the time we step to as `prevbefore` the time provided. +# Clipping at t1 saves us! We need to clip at at least 1 ULP. +def test_jump_at_t1_with_large_t1_in_float32(): + t0 = jnp.array(0.0, dtype=jnp.float32) + t1 = jnp.array(1e3, dtype=jnp.float32) + dt0 = jnp.array(0.01, dtype=jnp.float32) + y0 = jnp.array(1, dtype=jnp.float32) + saveat = diffrax.SaveAt(ts=t1[None]) + ssc = diffrax.ClipStepSizeController( + diffrax.PIDController(atol=1e-6, rtol=1e-6), jump_ts=t1[None] + ) + sol = diffrax.diffeqsolve( + diffrax.ODETerm(lambda t, y, args: -y), + diffrax.Heun(), + t0=t0, + t1=t1, + dt0=dt0, + y0=y0, + stepsize_controller=ssc, + saveat=saveat, + ) + assert sol.ts == jnp.array([t1]) From b7dc392f8d8748cdf3fcee9b962ca597de760be5 Mon Sep 17 00:00:00 2001 From: Owen Lockwood <42878312+lockwo@users.noreply.github.com> Date: Sun, 3 Aug 2025 17:54:38 -0400 Subject: [PATCH 13/24] adapt --- diffrax/_solver/align.py | 5 ++++- diffrax/_solver/spark.py | 4 ++-- 2 files changed, 6 insertions(+), 3 deletions(-) diff --git a/diffrax/_solver/align.py b/diffrax/_solver/align.py index 806a8a8f..750469b6 100644 --- a/diffrax/_solver/align.py +++ b/diffrax/_solver/align.py @@ -15,6 +15,7 @@ UnderdampedLangevinTuple, UnderdampedLangevinX, ) +from .base import AbstractAdaptiveSolver from .foster_langevin_srk import ( AbstractCoeffs, AbstractFosterLangevinSRK, @@ -44,7 +45,9 @@ def __init__(self, beta, a1, b1, aa, chh): _ErrorEstimate = UnderdampedLangevinTuple -class ALIGN(AbstractFosterLangevinSRK[_ALIGNCoeffs, _ErrorEstimate]): +class ALIGN( + AbstractFosterLangevinSRK[_ALIGNCoeffs, _ErrorEstimate], AbstractAdaptiveSolver +): r"""The Adaptive Langevin via Interpolated Gradients and Noise method designed by James Foster. This is a second order solver for the Underdamped Langevin Diffusion, and accepts terms of the form diff --git a/diffrax/_solver/spark.py b/diffrax/_solver/spark.py index bfaa0323..dda948b9 100644 --- a/diffrax/_solver/spark.py +++ b/diffrax/_solver/spark.py @@ -3,7 +3,7 @@ import equinox.internal as eqxi import numpy as np -from .base import AbstractStratonovichSolver +from .base import AbstractAdaptiveSolver, AbstractStratonovichSolver from .srk import AbstractSRK, GeneralCoeffs, StochasticButcherTableau @@ -35,7 +35,7 @@ ) -class SPaRK(AbstractSRK, AbstractStratonovichSolver): +class SPaRK(AbstractSRK, AbstractStratonovichSolver, AbstractAdaptiveSolver): r"""The Splitting Path Runge-Kutta method. It uses three evaluations of the drift and diffusion per step, and has the following From 8a72ee5cc662af11402100737b254e37002037dd Mon Sep 17 00:00:00 2001 From: Patrick Kidger <33688385+patrick-kidger@users.noreply.github.com> Date: Sun, 31 Aug 2025 13:47:19 +0200 Subject: [PATCH 14/24] Fixes 681 --- diffrax/_integrate.py | 10 ++-------- 1 file changed, 2 insertions(+), 8 deletions(-) diff --git a/diffrax/_integrate.py b/diffrax/_integrate.py index bc319d40..1855d4ae 100644 --- a/diffrax/_integrate.py +++ b/diffrax/_integrate.py @@ -161,10 +161,7 @@ def _check(term_cls, term, term_contr_kwargs, yi): pass elif n_term_args == 2: vf_type_expected, control_type_expected = term_args - try: - vf_type = eqx.filter_eval_shape(term.vf, t, yi, args) - except Exception as e: - raise ValueError(f"Error while tracing {term}.vf: " + str(e)) + vf_type = eqx.filter_eval_shape(term.vf, t, yi, args) vf_type_compatible = eqx.filter_eval_shape( better_isinstance, vf_type, vf_type_expected ) @@ -173,10 +170,7 @@ def _check(term_cls, term, term_contr_kwargs, yi): contr = ft.partial(term.contr, **term_contr_kwargs) # Work around https://github.com/google/jax/issues/21825 - try: - control_type = eqx.filter_eval_shape(contr, t, t) - except Exception as e: - raise ValueError(f"Error while tracing {term}.contr: " + str(e)) + control_type = eqx.filter_eval_shape(contr, t, t) control_type_compatible = eqx.filter_eval_shape( better_isinstance, control_type, control_type_expected ) From 2fd3ef34ef4267425b28b4c1338803dd1ec22e9a Mon Sep 17 00:00:00 2001 From: Patrick Kidger <33688385+patrick-kidger@users.noreply.github.com> Date: Fri, 12 Sep 2025 17:17:32 +0200 Subject: [PATCH 15/24] Added benchmarking FAQ --- docs/further_details/faq.md | 62 +++++++++++++++++++++++++++++++++---- 1 file changed, 56 insertions(+), 6 deletions(-) diff --git a/docs/further_details/faq.md b/docs/further_details/faq.md index fe30a402..79d44d6c 100644 --- a/docs/further_details/faq.md +++ b/docs/further_details/faq.md @@ -1,15 +1,65 @@ # FAQ -### Compilation is taking a long time. - -- Set `dt0=`, e.g. `diffeqsolve(..., dt0=0.01)`. In contrast `dt0=None` will determine the initial step size automatically, but will increase compilation time. -- Prefer `SaveAt(t0=True, t1=True)` over `SaveAt(ts=[t0, t1])`, if possible. -- It's an internal (subject-to-change) API, but you can also try adding `equinox.internal.noinline` to your vector field (s), e.g. `ODETerm(noinline(...))`. This stages the vector field out into a separate compilation graph. This can greatly decrease compilation time whilst greatly increasing runtime. - ### The solve is taking loads of steps / I'm getting NaN gradients / other weird behaviour. Try switching to 64-bit precision. (Instead of the 32-bit that is the default in JAX.) [See here](https://jax.readthedocs.io/en/latest/notebooks/Common_Gotchas_in_JAX.html#double-64bit-precision). +### Diffrax seem to be slower than ? + +Questions of this form are a fairly common source of issues in the Diffrax issue tracker! In practice, Diffrax is consistently amongst the fastest ODE solvers, and these usually stem from incorrect usage (e.g. recompiling your JAX program on each invocation) or comparisons (e.g. using different solvers/tolerances in each implementation). + +Here's a list of some of the things to keep in mind when performing such comparisons: + +1. First of all, the usual list of JAX profiling concerns: + + a. Make sure that your JAX program is compiled only once, and not repeatedly on each invocation (for example by passing in different raw Python floats each time). Use [`equinox.debug.assert_max_traces(max_traces=1)`](https://docs.kidger.site/equinox/api/debug/#equinox.debug.assert_max_traces) to debug this. + + b. Your entire computation should be wrapped in a single `jax.jit`'d function (or equivalently `equinox.filter_jit`). + + c. Run this function in advance (to JIT-compile it), before running it again to measure its speed. + + d. Make sure not to include any code that is ran outside of the JIT'd function in your timings. + + e. Make sure to call `jax.block_until_ready` on the output of the the function. + + Typically your code should follow this template: + ```python + import equinox as eqx + import jax + import timeit + + @jax.jit + @eqx.debug.assert_max_traces(max_traces=1) + def run(x): + ... + + x = ... + run(x) # compile + execution_time = min(timeit.repeat(lambda: jax.block_until_ready(run(x)), number=1, repeat=20)) + ``` + +2. Use the same ODE solver in both implementations to get an apples-to-apples comparison. It's not surprising that different solvers give different performance characteristics. (And if one implementation does not provide a solver that the other does, then no comparison can be made.) + +3. Use the same step size control in both implementations. + + a. If using adaptive step sizes then note that tolerances (the `rtol`, `atol` in `diffeqsolve(..., stepsize_controller=PIDController(rtol=..., atol=...))`) have solver- and implementation-specific meanings, so having these be the same is not enough. Aim to have roughly the same number of steps instead. You can check the number of steps taken in Diffrax via `diffeqsolve(...).stats['num_steps']`. + + b. If using an automatic initial step size (`diffeqsolve(..., dt0=None)`) then use this (or disable this) in both implementations. + +4. If comparing to other JAX implementations, then make sure to set `import os; os.environ["EQX_ON_ERROR"] = "nan"` at the top of your script (before you import Diffrax or Equinox). This will disable various runtime correctness checks performed by Diffrax that are are typically not performed by other JAX frameworks. These add a few milliseconds of overhead that typically does not matter in real-word usage but may be large enough to appear in microbenchmarks. + + a. If comparing to a loop-over-steps using `jax.lax.scan`, then the equivalent step size control in Diffrax is `diffeqsolve(..., stepsize_controller=StepTo(...))`. + +5. If you'd like to be really precise, then the best way to benchmark competing implementations is with a work-precision diagram: solve your ODE once with very tight tolerances and a very accurate solver (in any implementation). Then for each implementation: vary the tolerances or step sizes, and plot the time for the solve against and the numerical difference between the solution and the very accurate solution. This isn't required but is the gold-standard for benchmark comparisons. + +6. Both implementations should use the same precision (`float32` vs `float64`). Note that JAX defaults to 32-bit precision and requires a flag to enable 64-bit precision. + +7. The problem being solved should be large enough (ideally at least 100 milliseconds to solve) that you are not simply measuring various small overheads in different frameworks. + +Take a look at [Diffrax issue #82](https://github.com/patrick-kidger/diffrax/issues/82) for a good example of how seemingly-reasonable benchmarks can hide a few pitfalls! + +If you think you have a performance issue – after checking all of the above! – then feel free to open an issue on the Diffrax issue page. You should include a code snippet that demonstrates the issue; typically this should not be more than around 50 lines long if we are going to be able to volunteer to help you debug it :-). + ### How does this compare to `jax.experimental.ode.odeint`? The equivalent solver in Diffrax is: From 43f82dc5aa6db586e96349c057890ff8604bc333 Mon Sep 17 00:00:00 2001 From: Patrick Kidger <33688385+patrick-kidger@users.noreply.github.com> Date: Fri, 3 Oct 2025 14:21:36 +0200 Subject: [PATCH 16/24] Standardised infra --- .github/workflows/release.yml | 6 +++--- .github/workflows/run_tests.yml | 11 +++++++---- .gitignore | 1 + .pre-commit-config.yaml | 8 ++++---- CONTRIBUTING.md | 18 +++++------------- diffrax/_adjoint.py | 3 +-- diffrax/_brownian/path.py | 6 +++--- diffrax/_brownian/tree.py | 6 +++--- diffrax/_integrate.py | 6 +++++- diffrax/_solution.py | 10 +++++----- diffrax/_solver/dopri5.py | 6 +++--- diffrax/_solver/dopri8.py | 6 +++--- diffrax/_solver/euler.py | 6 +++--- diffrax/_solver/euler_heun.py | 6 +++--- diffrax/_solver/foster_langevin_srk.py | 6 +++--- diffrax/_solver/implicit_euler.py | 6 +++--- diffrax/_solver/kencarp3.py | 6 +++--- diffrax/_solver/kencarp4.py | 6 +++--- diffrax/_solver/kencarp5.py | 6 +++--- diffrax/_solver/leapfrog_midpoint.py | 6 +++--- diffrax/_solver/milstein.py | 12 ++++++------ diffrax/_solver/reversible_heun.py | 6 +++--- diffrax/_solver/runge_kutta.py | 7 ++++++- diffrax/_solver/semi_implicit_euler.py | 10 +++++----- diffrax/_solver/srk.py | 3 +-- diffrax/_solver/tsit5.py | 6 +++--- diffrax/_step_size_controller/clip.py | 3 +-- diffrax/_step_size_controller/pid.py | 2 +- diffrax/_term.py | 8 ++++---- diffrax/_typing.py | 2 +- docs/devdocs/srk_example.ipynb | 10 +++++----- pyproject.toml | 20 +++++++++++++++++--- test/helpers.py | 4 ++-- test/requirements.txt | 6 ------ test/test_brownian.py | 1 + test/test_saveat_solution.py | 4 ++++ test/test_solver.py | 6 +++--- test/test_underdamped_langevin.py | 12 ++++++------ 38 files changed, 136 insertions(+), 121 deletions(-) delete mode 100644 test/requirements.txt diff --git a/.github/workflows/release.yml b/.github/workflows/release.yml index 861187e0..82cc5c2f 100644 --- a/.github/workflows/release.yml +++ b/.github/workflows/release.yml @@ -10,14 +10,14 @@ jobs: runs-on: ubuntu-latest steps: - name: Release - uses: patrick-kidger/action_update_python_project@v6 + uses: patrick-kidger/action_update_python_project@v8 with: python-version: "3.11" test-script: | cp -r ${{ github.workspace }}/test ./test cp ${{ github.workspace }}/pyproject.toml ./pyproject.toml - python -m pip install -r ./test/requirements.txt - python -m test + uv sync --extra tests --no-install-project --inexact + uv run --no-sync pytest pypi-token: ${{ secrets.pypi_token }} github-user: patrick-kidger github-token: ${{ github.token }} diff --git a/.github/workflows/run_tests.yml b/.github/workflows/run_tests.yml index b209bb3d..0137d757 100644 --- a/.github/workflows/run_tests.yml +++ b/.github/workflows/run_tests.yml @@ -23,13 +23,16 @@ jobs: - name: Install dependencies run: | python -m pip install --upgrade pip - python -m pip install -r ./test/requirements.txt - + python -m pip install '.[dev,docs,tests]' - name: Checks with pre-commit - uses: pre-commit/action@v3.0.1 + run: | + pre-commit run --all-files - name: Test with pytest run: | - python -m pip install . python -m test + + - name: Check that documentation can be built. + run: | + mkdocs build diff --git a/.gitignore b/.gitignore index daf54d8c..176aa2af 100644 --- a/.gitignore +++ b/.gitignore @@ -8,3 +8,4 @@ site/ .pymon .idea/ .venv/ +uv.lock diff --git a/.pre-commit-config.yaml b/.pre-commit-config.yaml index 7f8df8ce..29b072c4 100644 --- a/.pre-commit-config.yaml +++ b/.pre-commit-config.yaml @@ -8,15 +8,15 @@ repos: files: ^pyproject\.toml$ additional_dependencies: ["toml-sort==0.23.1"] - repo: https://github.com/astral-sh/ruff-pre-commit - rev: v0.2.2 + rev: v0.13.0 hooks: - id: ruff-format # formatter - types_or: [ python, pyi, jupyter ] + types_or: [ python, pyi, jupyter, toml ] - id: ruff # linter - types_or: [ python, pyi, jupyter ] + types_or: [ python, pyi, jupyter, toml ] args: [ --fix ] - repo: https://github.com/RobertCraigie/pyright-python - rev: v1.1.350 + rev: v1.1.405 hooks: - id: pyright additional_dependencies: [equinox, jax, jaxtyping, optax, optimistix, lineax, pytest, typeguard==2.13.3, typing_extensions, wadler_lindig] diff --git a/CONTRIBUTING.md b/CONTRIBUTING.md index 1c9b3ced..78f188a3 100644 --- a/CONTRIBUTING.md +++ b/CONTRIBUTING.md @@ -8,23 +8,15 @@ Contributions (pull requests) are very welcome! Here's how to get started. First fork the library on GitHub. -Then clone and install the library in development mode: +Then clone and install the library: ```bash git clone https://github.com/your-username-here/diffrax.git cd diffrax -pip install -e . +pip install -e '.[dev]' +pre-commit install # `pre-commit` is installed by `pip` on the previous line ``` -Then install the pre-commit hook: - -```bash -pip install pre-commit -pre-commit install -``` - -These hooks use ruff to lint and format the code, and pyright to type-check it. - --- **If you're making changes to the code:** @@ -34,8 +26,8 @@ Now make your changes. Make sure to include additional tests if necessary. Next verify the tests all pass: ```bash -pip install -r test/requirements.txt -pytest +pip install -e '.[tests]' +pytest # `pytest` is installed by `pip` on the previous line. ``` Then push your changes back to your fork of the repository: diff --git a/diffrax/_adjoint.py b/diffrax/_adjoint.py index 69a59b7a..7bc081b9 100644 --- a/diffrax/_adjoint.py +++ b/diffrax/_adjoint.py @@ -362,8 +362,7 @@ def loop( if is_unsafe_sde(terms): kind = "lax" msg = ( - "Cannot reverse-mode autodifferentiate when using " - "`UnsafeBrownianPath`." + "Cannot reverse-mode autodifferentiate when using `UnsafeBrownianPath`." ) elif max_steps is None: kind = "lax" diff --git a/diffrax/_brownian/path.py b/diffrax/_brownian/path.py index f97eebf1..61f49644 100644 --- a/diffrax/_brownian/path.py +++ b/diffrax/_brownian/path.py @@ -62,9 +62,9 @@ class UnsafeBrownianPath(AbstractBrownianPath): """ shape: PyTree[jax.ShapeDtypeStruct] = eqx.field(static=True) - levy_area: type[ - BrownianIncrement | SpaceTimeLevyArea | SpaceTimeTimeLevyArea - ] = eqx.field(static=True) + levy_area: type[BrownianIncrement | SpaceTimeLevyArea | SpaceTimeTimeLevyArea] = ( + eqx.field(static=True) + ) key: PRNGKeyArray def __init__( diff --git a/diffrax/_brownian/tree.py b/diffrax/_brownian/tree.py index fd0ede84..8a430668 100644 --- a/diffrax/_brownian/tree.py +++ b/diffrax/_brownian/tree.py @@ -235,9 +235,9 @@ class VirtualBrownianTree(AbstractBrownianPath): t1: RealScalarLike tol: RealScalarLike shape: PyTree[jax.ShapeDtypeStruct] = eqx.field(static=True) - levy_area: type[ - BrownianIncrement | SpaceTimeLevyArea | SpaceTimeTimeLevyArea - ] = eqx.field(static=True) + levy_area: type[BrownianIncrement | SpaceTimeLevyArea | SpaceTimeTimeLevyArea] = ( + eqx.field(static=True) + ) key: PyTree[PRNGKeyArray] _spline: _Spline = eqx.field(static=True) diff --git a/diffrax/_integrate.py b/diffrax/_integrate.py index 1855d4ae..6fc38ce3 100644 --- a/diffrax/_integrate.py +++ b/diffrax/_integrate.py @@ -3,6 +3,7 @@ from collections.abc import Callable from typing import ( # noqa: UP035 Any, + cast, get_args, get_origin, Tuple, @@ -1164,7 +1165,10 @@ def _wrap(term): def _get_tols(x): outs = [] for attr in ("rtol", "atol", "norm"): - if getattr(solver.root_finder, attr) is use_stepsize_tol: + if ( + getattr(cast(AbstractImplicitSolver, solver).root_finder, attr) + is use_stepsize_tol + ): outs.append(getattr(x, attr)) return tuple(outs) diff --git a/diffrax/_solution.py b/diffrax/_solution.py index 3abe2725..392c447e 100644 --- a/diffrax/_solution.py +++ b/diffrax/_solution.py @@ -10,7 +10,7 @@ from ._path import AbstractPath -class RESULTS(optx.RESULTS): # pyright: ignore +class RESULTS(optx.RESULTS): # pyright: ignore[reportGeneralTypeIssues] successful = "" max_steps_reached = ( "The maximum number of solver steps was reached. Try increasing `max_steps`." @@ -121,8 +121,8 @@ class Solution(AbstractPath): # the structure of `subs`. # SaveAt(fn=...) means that `ys` will then follow with arbitrary sub-dependent # PyTree structures. - ts: PyTree[Real[Array, " ?times"], " S"] | None - ys: PyTree[Shaped[Array, "?times ?*shape"], "S ..."] | None + ts: PyTree[Real[Array, " ?times"], " S"] | None # pyright: ignore[reportUndefinedVariable] + ys: PyTree[Shaped[Array, "?times ?*shape"], "S ..."] | None # pyright: ignore interpolation: DenseInterpolation | None stats: dict[str, Any] result: RESULTS @@ -133,7 +133,7 @@ class Solution(AbstractPath): def evaluate( self, t0: RealScalarLike, t1: RealScalarLike | None = None, left: bool = True - ) -> PyTree[Shaped[Array, "?*shape"], " Y"]: + ) -> PyTree[Shaped[Array, "?*shape"], " Y"]: # pyright: ignore[reportUndefinedVariable] """If dense output was saved, then evaluate the solution at any point in the region of integration `self.t0` to `self.t1`. @@ -153,7 +153,7 @@ def evaluate( def derivative( self, t: RealScalarLike, left: bool = True - ) -> PyTree[Shaped[Array, "?*shape"], " Y"]: + ) -> PyTree[Shaped[Array, "?*shape"], " Y"]: # pyright: ignore[reportUndefinedVariable] r"""If dense output was saved, then calculate an **approximation** to the derivative of the solution at any point in the region of integration `self.t0` to `self.t1`. diff --git a/diffrax/_solver/dopri5.py b/diffrax/_solver/dopri5.py index 4a3cedfe..325f717f 100644 --- a/diffrax/_solver/dopri5.py +++ b/diffrax/_solver/dopri5.py @@ -91,9 +91,9 @@ class Dopri5(AbstractERK): """ tableau: ClassVar[ButcherTableau] = _dopri5_tableau - interpolation_cls: ClassVar[ - Callable[..., _Dopri5Interpolation] - ] = _Dopri5Interpolation + interpolation_cls: ClassVar[Callable[..., _Dopri5Interpolation]] = ( + _Dopri5Interpolation + ) def order(self, terms): del terms diff --git a/diffrax/_solver/dopri8.py b/diffrax/_solver/dopri8.py index 4801eccc..958d8819 100644 --- a/diffrax/_solver/dopri8.py +++ b/diffrax/_solver/dopri8.py @@ -340,9 +340,9 @@ class Dopri8(AbstractERK): """ tableau: ClassVar[ButcherTableau] = _dopri8_tableau - interpolation_cls: ClassVar[ - Callable[..., _Dopri8Interpolation] - ] = _Dopri8Interpolation + interpolation_cls: ClassVar[Callable[..., _Dopri8Interpolation]] = ( + _Dopri8Interpolation + ) def order(self, terms): del terms diff --git a/diffrax/_solver/euler.py b/diffrax/_solver/euler.py index 7ed11381..b1a323f7 100644 --- a/diffrax/_solver/euler.py +++ b/diffrax/_solver/euler.py @@ -24,9 +24,9 @@ class Euler(AbstractItoSolver): """ term_structure: ClassVar = AbstractTerm - interpolation_cls: ClassVar[ - Callable[..., LocalLinearInterpolation] - ] = LocalLinearInterpolation + interpolation_cls: ClassVar[Callable[..., LocalLinearInterpolation]] = ( + LocalLinearInterpolation + ) def order(self, terms): return 1 diff --git a/diffrax/_solver/euler_heun.py b/diffrax/_solver/euler_heun.py index 4940cfa5..97fbaee2 100644 --- a/diffrax/_solver/euler_heun.py +++ b/diffrax/_solver/euler_heun.py @@ -29,9 +29,9 @@ class EulerHeun(AbstractStratonovichSolver): term_structure: ClassVar = MultiTerm[ tuple[AbstractTerm[Any, RealScalarLike], AbstractTerm] ] - interpolation_cls: ClassVar[ - Callable[..., LocalLinearInterpolation] - ] = LocalLinearInterpolation + interpolation_cls: ClassVar[Callable[..., LocalLinearInterpolation]] = ( + LocalLinearInterpolation + ) def order(self, terms): return 1 diff --git a/diffrax/_solver/foster_langevin_srk.py b/diffrax/_solver/foster_langevin_srk.py index 717477d5..ea1026fa 100644 --- a/diffrax/_solver/foster_langevin_srk.py +++ b/diffrax/_solver/foster_langevin_srk.py @@ -231,9 +231,9 @@ def _choose(tay_leaf, direct_leaf): if inner is sentinel: inner = jtu.tree_structure(out) else: - assert ( - jtu.tree_structure(out) == inner - ), f"Expected {inner}, got {jtu.tree_structure(out)}" + assert jtu.tree_structure(out) == inner, ( + f"Expected {inner}, got {jtu.tree_structure(out)}" + ) return out diff --git a/diffrax/_solver/implicit_euler.py b/diffrax/_solver/implicit_euler.py index 064209da..68477c78 100644 --- a/diffrax/_solver/implicit_euler.py +++ b/diffrax/_solver/implicit_euler.py @@ -35,9 +35,9 @@ class ImplicitEuler(AbstractImplicitSolver, AbstractAdaptiveSolver): # # We don't use it as this seems to be quite a bad choice for low-order solvers: it # produces very oscillatory interpolations. - interpolation_cls: ClassVar[ - Callable[..., LocalLinearInterpolation] - ] = LocalLinearInterpolation + interpolation_cls: ClassVar[Callable[..., LocalLinearInterpolation]] = ( + LocalLinearInterpolation + ) root_finder: optx.AbstractRootFinder = with_stepsize_controller_tols(optx.Chord)() root_find_max_steps: int = 10 diff --git a/diffrax/_solver/kencarp3.py b/diffrax/_solver/kencarp3.py index f15c7c03..cda04489 100644 --- a/diffrax/_solver/kencarp3.py +++ b/diffrax/_solver/kencarp3.py @@ -163,9 +163,9 @@ class KenCarp3(AbstractRungeKutta, AbstractImplicitSolver): _explicit_tableau, _implicit_tableau ) calculate_jacobian: ClassVar[CalculateJacobian] = CalculateJacobian.second_stage - interpolation_cls: ClassVar[ - Callable[..., _KenCarp3Interpolation] - ] = _KenCarp3Interpolation + interpolation_cls: ClassVar[Callable[..., _KenCarp3Interpolation]] = ( + _KenCarp3Interpolation + ) root_finder: optx.AbstractRootFinder = with_stepsize_controller_tols(VeryChord)() root_find_max_steps: int = 10 diff --git a/diffrax/_solver/kencarp4.py b/diffrax/_solver/kencarp4.py index 4764900f..752cf38c 100644 --- a/diffrax/_solver/kencarp4.py +++ b/diffrax/_solver/kencarp4.py @@ -166,9 +166,9 @@ class KenCarp4(AbstractRungeKutta, AbstractImplicitSolver): _explicit_tableau, _implicit_tableau ) calculate_jacobian: ClassVar[CalculateJacobian] = CalculateJacobian.second_stage - interpolation_cls: ClassVar[ - Callable[..., _KenCarp4Interpolation] - ] = _KenCarp4Interpolation + interpolation_cls: ClassVar[Callable[..., _KenCarp4Interpolation]] = ( + _KenCarp4Interpolation + ) root_finder: optx.AbstractRootFinder = with_stepsize_controller_tols(VeryChord)() root_find_max_steps: int = 10 diff --git a/diffrax/_solver/kencarp5.py b/diffrax/_solver/kencarp5.py index ba9af78c..b5b0f213 100644 --- a/diffrax/_solver/kencarp5.py +++ b/diffrax/_solver/kencarp5.py @@ -233,9 +233,9 @@ class KenCarp5(AbstractRungeKutta, AbstractImplicitSolver): _explicit_tableau, _implicit_tableau ) calculate_jacobian: ClassVar[CalculateJacobian] = CalculateJacobian.second_stage - interpolation_cls: ClassVar[ - Callable[..., _KenCarp5Interpolation] - ] = _KenCarp5Interpolation + interpolation_cls: ClassVar[Callable[..., _KenCarp5Interpolation]] = ( + _KenCarp5Interpolation + ) root_finder: optx.AbstractRootFinder = with_stepsize_controller_tols(VeryChord)() root_find_max_steps: int = 10 diff --git a/diffrax/_solver/leapfrog_midpoint.py b/diffrax/_solver/leapfrog_midpoint.py index 76487dfd..ddcaa12e 100644 --- a/diffrax/_solver/leapfrog_midpoint.py +++ b/diffrax/_solver/leapfrog_midpoint.py @@ -44,9 +44,9 @@ class LeapfrogMidpoint(AbstractSolver): """ term_structure: ClassVar = AbstractTerm - interpolation_cls: ClassVar[ - Callable[..., LocalLinearInterpolation] - ] = LocalLinearInterpolation + interpolation_cls: ClassVar[Callable[..., LocalLinearInterpolation]] = ( + LocalLinearInterpolation + ) def order(self, terms): return 2 diff --git a/diffrax/_solver/milstein.py b/diffrax/_solver/milstein.py index 69893716..897bc9c1 100644 --- a/diffrax/_solver/milstein.py +++ b/diffrax/_solver/milstein.py @@ -44,9 +44,9 @@ class StratonovichMilstein(AbstractStratonovichSolver): term_structure: ClassVar = MultiTerm[ tuple[AbstractTerm[Any, RealScalarLike], AbstractTerm] ] - interpolation_cls: ClassVar[ - Callable[..., LocalLinearInterpolation] - ] = LocalLinearInterpolation + interpolation_cls: ClassVar[Callable[..., LocalLinearInterpolation]] = ( + LocalLinearInterpolation + ) def order(self, terms): raise ValueError("`StratonovichMilstein` should not be used to solve ODEs.") @@ -123,9 +123,9 @@ class ItoMilstein(AbstractItoSolver): term_structure: ClassVar = MultiTerm[ tuple[AbstractTerm[Any, RealScalarLike], AbstractTerm] ] - interpolation_cls: ClassVar[ - Callable[..., LocalLinearInterpolation] - ] = LocalLinearInterpolation + interpolation_cls: ClassVar[Callable[..., LocalLinearInterpolation]] = ( + LocalLinearInterpolation + ) def order(self, terms): raise ValueError("`ItoMilstein` should not be used to solve ODEs.") diff --git a/diffrax/_solver/reversible_heun.py b/diffrax/_solver/reversible_heun.py index 9288d00d..91617d4f 100644 --- a/diffrax/_solver/reversible_heun.py +++ b/diffrax/_solver/reversible_heun.py @@ -36,9 +36,9 @@ class ReversibleHeun(AbstractAdaptiveSolver, AbstractStratonovichSolver): """ term_structure: ClassVar = AbstractTerm - interpolation_cls: ClassVar[ - Callable[..., LocalLinearInterpolation] - ] = LocalLinearInterpolation # TODO use something better than this? + interpolation_cls: ClassVar[Callable[..., LocalLinearInterpolation]] = ( + LocalLinearInterpolation # TODO use something better than this? + ) def order(self, terms): return 2 diff --git a/diffrax/_solver/runge_kutta.py b/diffrax/_solver/runge_kutta.py index f46d2045..9473ab44 100644 --- a/diffrax/_solver/runge_kutta.py +++ b/diffrax/_solver/runge_kutta.py @@ -358,6 +358,11 @@ class AbstractRungeKutta(AbstractAdaptiveSolver[_SolverState]): tableau: AbstractClassVar[ButcherTableau | MultiButcherTableau] calculate_jacobian: AbstractClassVar[CalculateJacobian] + if TYPE_CHECKING: + # Pretend that we're implicit + root_finder: ClassVar[optx.AbstractRootFinder] + root_find_max_steps: ClassVar[int] + def __init_subclass__(cls, **kwargs): super().__init_subclass__(**kwargs) if hasattr(cls, "tableau"): # Abstract subclasses may not have a tableau @@ -804,7 +809,7 @@ def embed_c(tab): ) implicit_predictor = np.zeros( (num_stages, num_stages), - dtype=np.result_type(*implicit_tableau.a_predictor), + dtype=np.result_type(*cast(tuple, implicit_tableau.a_predictor)), ) for i, a_predictor_i in enumerate(implicit_tableau.a_predictor): # pyright: ignore implicit_predictor[i + 1, : i + 1] = a_predictor_i diff --git a/diffrax/_solver/semi_implicit_euler.py b/diffrax/_solver/semi_implicit_euler.py index 9e8a92ed..34122fe3 100644 --- a/diffrax/_solver/semi_implicit_euler.py +++ b/diffrax/_solver/semi_implicit_euler.py @@ -14,8 +14,8 @@ _ErrorEstimate: TypeAlias = None _SolverState: TypeAlias = None -Ya: TypeAlias = PyTree[Float[ArrayLike, "?*y"], " Y"] -Yb: TypeAlias = PyTree[Float[ArrayLike, "?*y"], " Y"] +Ya: TypeAlias = PyTree[Float[ArrayLike, "?*y"], " Y"] # pyright: ignore[reportUndefinedVariable] +Yb: TypeAlias = PyTree[Float[ArrayLike, "?*y"], " Y"] # pyright: ignore[reportUndefinedVariable] class SemiImplicitEuler(AbstractSolver): @@ -26,9 +26,9 @@ class SemiImplicitEuler(AbstractSolver): """ term_structure: ClassVar = (AbstractTerm, AbstractTerm) - interpolation_cls: ClassVar[ - Callable[..., LocalLinearInterpolation] - ] = LocalLinearInterpolation + interpolation_cls: ClassVar[Callable[..., LocalLinearInterpolation]] = ( + LocalLinearInterpolation + ) def order(self, terms): return 1 diff --git a/diffrax/_solver/srk.py b/diffrax/_solver/srk.py index 1af9579f..ddb03095 100644 --- a/diffrax/_solver/srk.py +++ b/diffrax/_solver/srk.py @@ -54,8 +54,7 @@ class AbstractStochasticCoeffs(eqx.Module): b_error: eqx.AbstractVar[Float[np.ndarray, " s"] | None] @abc.abstractmethod - def check(self) -> int: - ... + def check(self) -> int: ... class AdditiveCoeffs(AbstractStochasticCoeffs): diff --git a/diffrax/_solver/tsit5.py b/diffrax/_solver/tsit5.py index 3060088a..7dc7a14f 100644 --- a/diffrax/_solver/tsit5.py +++ b/diffrax/_solver/tsit5.py @@ -181,9 +181,9 @@ class Tsit5(AbstractERK): """ tableau: ClassVar[ButcherTableau] = _tsit5_tableau - interpolation_cls: ClassVar[ - Callable[..., _Tsit5Interpolation] - ] = _Tsit5Interpolation + interpolation_cls: ClassVar[Callable[..., _Tsit5Interpolation]] = ( + _Tsit5Interpolation + ) def order(self, terms): return 5 diff --git a/diffrax/_step_size_controller/clip.py b/diffrax/_step_size_controller/clip.py index c8a5cb76..0a642d6e 100644 --- a/diffrax/_step_size_controller/clip.py +++ b/diffrax/_step_size_controller/clip.py @@ -220,8 +220,7 @@ def __init__( self.jump_ts = _none_or_sorted_array(jump_ts) if (store_rejected_steps is not None) and (store_rejected_steps <= 0): raise ValueError( - "`store_rejected_steps must either be `None`" - " or a non-negative integer." + "`store_rejected_steps must either be `None` or a non-negative integer." ) self.store_rejected_steps = store_rejected_steps self.callback_on_reject = _callback_on_reject diff --git a/diffrax/_step_size_controller/pid.py b/diffrax/_step_size_controller/pid.py index 7fb034fb..1092f184 100644 --- a/diffrax/_step_size_controller/pid.py +++ b/diffrax/_step_size_controller/pid.py @@ -89,7 +89,7 @@ def intermediate(carry): # PIDController(... step_ts=s, jump_ts=j) this should return a # ClipStepSizeController(PIDController(...), s, j). class _MetaPID(type(eqx.Module)): - def __call__(cls, *args, **kwargs): # pyright: ignore[reportSelfClsParameterName] + def __call__(cls, *args, **kwargs): step_ts = kwargs.pop("step_ts", None) jump_ts = kwargs.pop("jump_ts", None) if step_ts is not None or jump_ts is not None: diff --git a/diffrax/_term.py b/diffrax/_term.py index 41f7af09..9b0c0314 100644 --- a/diffrax/_term.py +++ b/diffrax/_term.py @@ -3,7 +3,7 @@ import typing import warnings from collections.abc import Callable -from typing import cast, Generic, TypeAlias, TypeVar +from typing import Any, cast, Generic, TypeAlias, TypeVar import equinox as eqx import jax @@ -12,7 +12,7 @@ import lineax as lx import numpy as np from equinox.internal import ω -from jaxtyping import Array, ArrayLike, PyTree, PyTreeDef, Shaped +from jaxtyping import Array, ArrayLike, PyTree, Shaped from ._brownian import AbstractBrownianPath from ._custom_types import ( @@ -835,7 +835,7 @@ def _fn(_control): jac = make_jac(_fn)(control) assert vf_prod_tree is not sentinel - vf_prod_tree = cast(PyTreeDef, vf_prod_tree) + vf_prod_tree = cast(Any, vf_prod_tree) if jtu.tree_structure(None) in (vf_prod_tree, control_tree): # An unusual/not-useful edge case to handle. raise NotImplementedError( @@ -868,7 +868,7 @@ def _get_vf_tree(_, tree): jtu.tree_map(_get_vf_tree, control, vf) assert vf_prod_tree is not sentinel - vf_prod_tree = cast(PyTreeDef, vf_prod_tree) + vf_prod_tree = cast(Any, vf_prod_tree) vf = jtu.tree_transpose(control_tree, vf_prod_tree, vf) diff --git a/diffrax/_typing.py b/diffrax/_typing.py index 74627847..90a0b7b8 100644 --- a/diffrax/_typing.py +++ b/diffrax/_typing.py @@ -184,7 +184,7 @@ def _get_args_of_impl( if len(params) == 0: error_cls = cls else: - error_cls = cls[params] + error_cls = cls[params] # pyright: ignore[reportIndexIssue] raise TypeError( f"{error_cls} inherits from {base_cls} in multiple incompatible ways." ) diff --git a/docs/devdocs/srk_example.ipynb b/docs/devdocs/srk_example.ipynb index 39364def..319beee2 100644 --- a/docs/devdocs/srk_example.ipynb +++ b/docs/devdocs/srk_example.ipynb @@ -55,11 +55,6 @@ "source": [ "%env JAX_PLATFORM_NAME=cuda\n", "\n", - "from test.helpers import (\n", - " get_mlp_sde,\n", - " get_time_sde,\n", - " simple_sde_order,\n", - ")\n", "from warnings import simplefilter\n", "\n", "import diffrax\n", @@ -76,6 +71,11 @@ " SRA1,\n", ")\n", "from jax import config\n", + "from test.helpers import (\n", + " get_mlp_sde,\n", + " get_time_sde,\n", + " simple_sde_order,\n", + ")\n", "\n", "\n", "simplefilter(\"ignore\", category=FutureWarning)\n", diff --git a/pyproject.toml b/pyproject.toml index e7326e29..c7ec6d27 100644 --- a/pyproject.toml +++ b/pyproject.toml @@ -25,13 +25,15 @@ keywords = ["jax", "dynamical-systems", "differential-equations", "deep-learning license = {file = "LICENSE"} name = "diffrax" readme = "README.md" -requires-python = ">=3.10,<4.0" +requires-python = ">=3.10" urls = {repository = "https://github.com/patrick-kidger/diffrax"} version = "0.7.0" [project.optional-dependencies] +dev = ["pre-commit"] docs = [ "hippogriffe==0.2.2", + "griffe==1.7.3", "mkdocs==1.6.1", "mkdocs-include-exclude-files==0.1.0", "mkdocs-ipynb==0.1.1", @@ -40,6 +42,14 @@ docs = [ "mkdocstrings-python==1.16.8", "pymdown-extensions==10.14.3" ] +tests = [ + "beartype", + "jaxlib", + "optax", + "pytest", + "scipy", + "tqdm" +] [tool.hatch.build] include = ["diffrax/*"] @@ -60,10 +70,14 @@ src = [] [tool.ruff.lint] fixable = ["I001", "F401", "UP"] -ignore = ["E402", "E721", "E731", "E741", "F722", "UP038"] -ignore-init-module-imports = true +ignore = ["E402", "E721", "E731", "E741", "F722"] select = ["E", "F", "I001", "UP"] +[tool.ruff.lint.flake8-import-conventions.extend-aliases] +"collections" = "co" +"functools" = "ft" +"itertools" = "it" + [tool.ruff.lint.isort] combine-as-imports = true extra-standard-library = ["typing_extensions"] diff --git a/test/helpers.py b/test/helpers.py index b6311065..97b0f074 100644 --- a/test/helpers.py +++ b/test/helpers.py @@ -95,8 +95,8 @@ def tree_allclose(x, y, *, rtol=1e-5, atol=1e-8, equal_nan=False): def path_l2_dist( - ys1: PyTree[Shaped[Array, "repeats times ?*channels"], " T"], - ys2: PyTree[Shaped[Array, "repeats times ?*channels"], " T"], + ys1: PyTree[Shaped[Array, "repeats times ?*channels"], " T"], # pyright: ignore[reportUndefinedVariable] + ys2: PyTree[Shaped[Array, "repeats times ?*channels"], " T"], # pyright: ignore[reportUndefinedVariable] ): # first compute the square of the difference and sum over # all but the first two axes (which represent the number of samples diff --git a/test/requirements.txt b/test/requirements.txt deleted file mode 100644 index 9de88eb6..00000000 --- a/test/requirements.txt +++ /dev/null @@ -1,6 +0,0 @@ -beartype -jaxlib -optax -pytest -scipy -tqdm diff --git a/test/test_brownian.py b/test/test_brownian.py index 361c761d..d33bbda8 100644 --- a/test/test_brownian.py +++ b/test/test_brownian.py @@ -173,6 +173,7 @@ def _eval(key): else: w = values + assert isinstance(w, jax.Array) assert w.shape == (num_samples,) ref_dist = stats.norm(loc=0, scale=math.sqrt(dt)) _, pval = stats.kstest(w, ref_dist.cdf) diff --git a/test/test_saveat_solution.py b/test/test_saveat_solution.py index 3db4e7ba..7b8b9dfc 100644 --- a/test/test_saveat_solution.py +++ b/test/test_saveat_solution.py @@ -1,5 +1,6 @@ import contextlib import math +from typing import cast import diffrax import equinox as eqx @@ -119,6 +120,7 @@ def test_saveat_solution(): assert sol.ts.shape == (4096,) # pyright: ignore assert sol.ys.shape == (4096, 1) # pyright: ignore _ts = jnp.where(sol.ts == jnp.inf, jnp.nan, sol.ts) + _ts = cast(jax.Array, _ts) with jax.numpy_rank_promotion("allow"): _ys = _y0 * jnp.exp(-0.5 * (_ts - _t0))[:, None] _ys = jnp.where(jnp.isnan(_ys), jnp.inf, _ys) @@ -140,6 +142,7 @@ def test_saveat_solution(): assert sol.ts.shape == (n,) # pyright: ignore assert sol.ys.shape == (n, 1) # pyright: ignore _ts = jnp.where(sol.ts == jnp.inf, jnp.nan, sol.ts) + _ts = cast(jax.Array, _ts) with jax.numpy_rank_promotion("allow"): _ys = _y0 * jnp.exp(-0.5 * (_ts - _t0))[:, None] _ys = jnp.where(jnp.isnan(_ys), jnp.inf, _ys) @@ -161,6 +164,7 @@ def test_saveat_solution(): assert sol.ts.shape == (n,) # pyright: ignore assert sol.ys.shape == (n, 1) # pyright: ignore _ts = jnp.where(sol.ts == jnp.inf, jnp.nan, sol.ts) + _ts = cast(jax.Array, _ts) with jax.numpy_rank_promotion("allow"): _ys = _y0 * jnp.exp(-0.5 * (_ts - _t0))[:, None] _ys = jnp.where(jnp.isnan(_ys), jnp.inf, _ys) diff --git a/test/test_solver.py b/test/test_solver.py index aa618712..a022f644 100644 --- a/test/test_solver.py +++ b/test/test_solver.py @@ -58,9 +58,9 @@ class _DoubleDopri5(diffrax.AbstractRungeKutta): tableau: ClassVar[diffrax.MultiButcherTableau] = diffrax.MultiButcherTableau( diffrax.Dopri5.tableau, diffrax.Dopri5.tableau ) - calculate_jacobian: ClassVar[ - diffrax.CalculateJacobian - ] = diffrax.CalculateJacobian.never + calculate_jacobian: ClassVar[diffrax.CalculateJacobian] = ( + diffrax.CalculateJacobian.never + ) @staticmethod def interpolation_cls(**kwargs): diff --git a/test/test_underdamped_langevin.py b/test/test_underdamped_langevin.py index 246506bb..53140fba 100644 --- a/test/test_underdamped_langevin.py +++ b/test/test_underdamped_langevin.py @@ -100,9 +100,9 @@ def test_shape(solver, dtype): # check that the output has the correct pytree structure and shape def check_shape(y0_leaf, sol_leaf): - assert ( - sol_leaf.shape == (7,) + y0_leaf.shape - ), f"shape={sol_leaf.shape}, expected={(7,) + y0_leaf.shape}" + assert sol_leaf.shape == (7,) + y0_leaf.shape, ( + f"shape={sol_leaf.shape}, expected={(7,) + y0_leaf.shape}" + ) assert sol_leaf.dtype == dtype, f"dtype={sol_leaf.dtype}, expected={dtype}" jtu.tree_map(check_shape, sde.y0, sol.ys) @@ -193,9 +193,9 @@ def get_dt_and_controller(level): ref_solution=true_sol, ) - assert ( - -0.2 < order - theoretical_order < 0.25 - ), f"order={order}, theoretical_order={theoretical_order}" + assert -0.2 < order - theoretical_order < 0.25, ( + f"order={order}, theoretical_order={theoretical_order}" + ) @pytest.mark.parametrize("solver_cls", _only_uld_solvers_cls()) From 6694c864ee59924305cc08a5a1d84bc7de57ed10 Mon Sep 17 00:00:00 2001 From: Patrick Kidger <33688385+patrick-kidger@users.noreply.github.com> Date: Fri, 3 Oct 2025 19:29:02 +0200 Subject: [PATCH 17/24] Fix failing test --- test/test_sde1.py | 3 ++- 1 file changed, 2 insertions(+), 1 deletion(-) diff --git a/test/test_sde1.py b/test/test_sde1.py index b50d014f..ad7318e6 100644 --- a/test/test_sde1.py +++ b/test/test_sde1.py @@ -157,7 +157,8 @@ def test_sde_strong_limit( ts_coarse = jnp.linspace(t0, t1, 2**level_coarse + 1, endpoint=True) contr_fine = diffrax.StepTo(ts=ts_fine) contr_coarse = diffrax.StepTo(ts=ts_coarse) - save_ts = jnp.linspace(t0, t1, 2**5 + 1, endpoint=True) + save_ts = ts_coarse[:: 2 ** (level_coarse - 5)] + assert len(save_ts) == 2**5 + 1 assert len(jnp.intersect1d(ts_fine, save_ts)) == len(save_ts) assert len(jnp.intersect1d(ts_coarse, save_ts)) == len(save_ts) saveat = diffrax.SaveAt(ts=save_ts) From 02d6b8a0a1c838a6d3d109418eec19aaac07e57e Mon Sep 17 00:00:00 2001 From: Philip Wijesinghe Date: Wed, 5 Nov 2025 11:46:58 +0000 Subject: [PATCH 18/24] fix float error in prev_dt step calculation that led to an infinite loop When: dt is clipped to dtmin, and we wish to continue solver (force_dtmin=True) Calculating if a step should be kept from: prev_dt = t1 - t0 (next_t1 = next_t0 + dt (in previous step)) keep_step = keep_step | (prev_dt <= self.dtmin) can result in float error for high t0 where prev_dt is never <= self.dtmin, and further steps are never accepted -> infinite loop Fix: add a keep_next_step: bool flag to controller_state, and track when we are, and continue to be, at dtmin --- diffrax/_step_size_controller/pid.py | 15 ++++++++++----- 1 file changed, 10 insertions(+), 5 deletions(-) diff --git a/diffrax/_step_size_controller/pid.py b/diffrax/_step_size_controller/pid.py index 1092f184..56c19f61 100644 --- a/diffrax/_step_size_controller/pid.py +++ b/diffrax/_step_size_controller/pid.py @@ -81,8 +81,8 @@ def intermediate(carry): return jnp.minimum(100 * h0, h1) -# _PidState = (prev_inv_scaled_error, prev_prev_inv_scaled_error) -_PidState = tuple[RealScalarLike, RealScalarLike] +# _PidState = (prev_inv_scaled_error, prev_prev_inv_scaled_error, keep_next_step) +_PidState = tuple[RealScalarLike, RealScalarLike, BoolScalarLike] # We use a metaclass for backwards compatibility. When a user calls @@ -388,6 +388,7 @@ def init( return t1, ( jnp.array(1.0, dtype=real_dtype), jnp.array(1.0, dtype=real_dtype), + False, ) def adapt_step_size( @@ -469,6 +470,7 @@ def adapt_step_size( ( prev_inv_scaled_error, prev_prev_inv_scaled_error, + keep_next_step, ) = controller_state error_order = self._get_error_order(error_order) prev_dt = t1 - t0 @@ -489,9 +491,9 @@ def _scale(_y0, _y1_candidate, _y_error): scaled_error = self.norm(jtu.tree_map(_scale, y0, y1_candidate, y_error)) keep_step = scaled_error < 1 - # Automatically keep the step if we're at dtmin. + # Automatically keep the step if it was at dtmin. if self.dtmin is not None: - keep_step = keep_step | (prev_dt <= self.dtmin) + keep_step = keep_step | keep_next_step # Make sure it's not a Python scalar and thus getting a ZeroDivisionError. inv_scaled_error = 1 / jnp.asarray(scaled_error) inv_scaled_error = lax.stop_gradient( @@ -545,6 +547,9 @@ def _scale(_y0, _y1_candidate, _y_error): if self.dtmin is not None: if not self.force_dtmin: result = RESULTS.where(dt < self.dtmin, RESULTS.dt_min_reached, result) + # flag next step to be kept if dtmin is reached + # or if it was reached previously and dt is unchanged + keep_next_step = (dt <= self.dtmin) | (keep_next_step & (factor == 1)) dt = jnp.maximum(dt, self.dtmin) next_t0 = jnp.where(keep_step, t1, t0) @@ -554,7 +559,7 @@ def _scale(_y0, _y1_candidate, _y_error): prev_inv_scaled_error = jnp.where( keep_step, prev_inv_scaled_error, prev_prev_inv_scaled_error ) - controller_state = inv_scaled_error, prev_inv_scaled_error + controller_state = inv_scaled_error, prev_inv_scaled_error, keep_next_step # made_jump is handled by ClipStepSizeController, so we automatically set it to # False return keep_step, next_t0, next_t1, False, controller_state, result From b91138f5f54880537ff036a612da97b5357330f1 Mon Sep 17 00:00:00 2001 From: Philip Wijesinghe Date: Thu, 6 Nov 2025 09:08:57 +0000 Subject: [PATCH 19/24] avoids accumulation of float precision errors in dt this solution makes sure that dt is reset to the desired dtmin value if the previous step was at dtmin and dt is unchanged (factor=1) if we do not reset dt then the recalculation of prev_dt = t1 - t0 will keep accumulating float precision errors with potential to drift away from the desired dtmin until a step that warrant a relaxation of step size (factor>1) these errors are likely to be minor, but i believe this is the intended behaviour --- diffrax/_step_size_controller/pid.py | 16 +++++++++------- 1 file changed, 9 insertions(+), 7 deletions(-) diff --git a/diffrax/_step_size_controller/pid.py b/diffrax/_step_size_controller/pid.py index 56c19f61..c0ccfbec 100644 --- a/diffrax/_step_size_controller/pid.py +++ b/diffrax/_step_size_controller/pid.py @@ -81,7 +81,7 @@ def intermediate(carry): return jnp.minimum(100 * h0, h1) -# _PidState = (prev_inv_scaled_error, prev_prev_inv_scaled_error, keep_next_step) +# _PidState = (prev_inv_scaled_error, prev_prev_inv_scaled_error, at_dtmin) _PidState = tuple[RealScalarLike, RealScalarLike, BoolScalarLike] @@ -470,7 +470,7 @@ def adapt_step_size( ( prev_inv_scaled_error, prev_prev_inv_scaled_error, - keep_next_step, + at_dtmin, ) = controller_state error_order = self._get_error_order(error_order) prev_dt = t1 - t0 @@ -493,7 +493,7 @@ def _scale(_y0, _y1_candidate, _y_error): keep_step = scaled_error < 1 # Automatically keep the step if it was at dtmin. if self.dtmin is not None: - keep_step = keep_step | keep_next_step + keep_step = keep_step | at_dtmin # Make sure it's not a Python scalar and thus getting a ZeroDivisionError. inv_scaled_error = 1 / jnp.asarray(scaled_error) inv_scaled_error = lax.stop_gradient( @@ -547,9 +547,11 @@ def _scale(_y0, _y1_candidate, _y_error): if self.dtmin is not None: if not self.force_dtmin: result = RESULTS.where(dt < self.dtmin, RESULTS.dt_min_reached, result) - # flag next step to be kept if dtmin is reached - # or if it was reached previously and dt is unchanged - keep_next_step = (dt <= self.dtmin) | (keep_next_step & (factor == 1)) + # if we are already at dtmin and dt is unchanged (factor == 1), + # reset dt to dtmin to avoid accumulating float precision errors + dt = jnp.where(at_dtmin & (factor == 1), self.dtmin, dt) + # this flags the next loop to accept step + at_dtmin = dt <= self.dtmin dt = jnp.maximum(dt, self.dtmin) next_t0 = jnp.where(keep_step, t1, t0) @@ -559,7 +561,7 @@ def _scale(_y0, _y1_candidate, _y_error): prev_inv_scaled_error = jnp.where( keep_step, prev_inv_scaled_error, prev_prev_inv_scaled_error ) - controller_state = inv_scaled_error, prev_inv_scaled_error, keep_next_step + controller_state = inv_scaled_error, prev_inv_scaled_error, at_dtmin # made_jump is handled by ClipStepSizeController, so we automatically set it to # False return keep_step, next_t0, next_t1, False, controller_state, result From 62bf87692e3b1a6e9b34b2c4398424e44d2cc8c7 Mon Sep 17 00:00:00 2001 From: Patrick Kidger <33688385+patrick-kidger@users.noreply.github.com> Date: Mon, 22 Dec 2025 17:35:15 +0100 Subject: [PATCH 20/24] Fixed case in which t0 is prevbefore a jump time --- .github/workflows/run_tests.yml | 2 +- benchmarks/against_scan.py | 2 +- diffrax/_step_size_controller/clip.py | 20 ++++++++++++++++- diffrax/_step_size_controller/pid.py | 2 +- pyproject.toml | 2 +- test/test_adaptive_stepsize_controller.py | 27 +++++++++++++++++++++++ 6 files changed, 50 insertions(+), 5 deletions(-) diff --git a/.github/workflows/run_tests.yml b/.github/workflows/run_tests.yml index 0137d757..5268212b 100644 --- a/.github/workflows/run_tests.yml +++ b/.github/workflows/run_tests.yml @@ -7,7 +7,7 @@ jobs: run-tests: strategy: matrix: - python-version: [ "3.10", "3.12" ] + python-version: [ "3.11", "3.13" ] os: [ ubuntu-latest ] fail-fast: false runs-on: ${{ matrix.os }} diff --git a/benchmarks/against_scan.py b/benchmarks/against_scan.py index 6c67655f..a7771c4e 100644 --- a/benchmarks/against_scan.py +++ b/benchmarks/against_scan.py @@ -36,7 +36,7 @@ def speedtest(fn, name): # INTEGRATE WITH scan -@jax.checkpoint # pyright: ignore +@jax.checkpoint def body(carry, t): u, v, dt = carry u = u + du(t, v, None) * dt diff --git a/diffrax/_step_size_controller/clip.py b/diffrax/_step_size_controller/clip.py index 0a642d6e..cac6ec1a 100644 --- a/diffrax/_step_size_controller/clip.py +++ b/diffrax/_step_size_controller/clip.py @@ -356,7 +356,7 @@ def callback(_keep_step, _t1): step_info = None else: step_index, step_ts = controller_state.step_info - # We actaully bump `next_t0` past any `step_ts` whilst checking where to + # We actually bump `next_t0` past any `step_ts` whilst checking where to # clip `next_t1`. This is in case we have a set up like the following: # ```python # ClipStepSizeController( @@ -376,6 +376,24 @@ def callback(_keep_step, _t1): else: jump_index, jump_ts = controller_state.jump_info next_t0, made_jump2 = _bump_next_t0(next_t0, jump_ts) + # This next line is to fix + # https://github.com/patrick-kidger/diffrax/issues/713 + # TODO: should we add this to the `step_ts` branch as well? + # + # What's going on here is that we may have + # the `diffeqsolve(t0=...)` be prevbefore a jump time (for example due to a + # previous diffeqsolve targeting that time), in which case during `.init` + # we will obtain `t0 = t1 = prevbefore(jump_time)`. + # The `_bump_next_t0` will then move `next_t0` to after the `jump_time`... + # whilst leaving `next_t1` unchanged! We actually end up `next_t1 < next_t0` + # which is very not okay. + # + # The fix is to ensure that `next_t1` is itself bumped to at least this + # value. As a final detail, we need to make it `nextafter` so that we don't + # have a zero-length interval – in this case an underlying PID controller + # would just never change the interval size at all, since it acts + # multiplicatively. (And even just 1 ULP is enough to unstick it.) + next_t1 = jnp.maximum(eqxi.nextafter(next_t0), next_t1) made_jump = made_jump | made_jump2 jump_index = _find_idx_with_hint(next_t0, jump_ts, jump_index) next_t1 = _clip_t(next_t1, jump_index, jump_ts, True) diff --git a/diffrax/_step_size_controller/pid.py b/diffrax/_step_size_controller/pid.py index c0ccfbec..752b9bed 100644 --- a/diffrax/_step_size_controller/pid.py +++ b/diffrax/_step_size_controller/pid.py @@ -93,7 +93,7 @@ def __call__(cls, *args, **kwargs): step_ts = kwargs.pop("step_ts", None) jump_ts = kwargs.pop("jump_ts", None) if step_ts is not None or jump_ts is not None: - return ClipStepSizeController(cls(*args, **kwargs), step_ts, jump_ts) + return ClipStepSizeController(cls(*args, **kwargs), step_ts, jump_ts) # pyright: ignore return super().__call__(*args, **kwargs) diff --git a/pyproject.toml b/pyproject.toml index c7ec6d27..459deb3c 100644 --- a/pyproject.toml +++ b/pyproject.toml @@ -25,7 +25,7 @@ keywords = ["jax", "dynamical-systems", "differential-equations", "deep-learning license = {file = "LICENSE"} name = "diffrax" readme = "README.md" -requires-python = ">=3.10" +requires-python = ">=3.11" urls = {repository = "https://github.com/patrick-kidger/diffrax"} version = "0.7.0" diff --git a/test/test_adaptive_stepsize_controller.py b/test/test_adaptive_stepsize_controller.py index 21c24e4a..d785d16e 100644 --- a/test/test_adaptive_stepsize_controller.py +++ b/test/test_adaptive_stepsize_controller.py @@ -7,6 +7,7 @@ import jax.numpy as jnp import jax.random as jr import jax.tree_util as jtu +import optimistix as optx import pytest from diffrax._step_size_controller.clip import _find_idx_with_hint from jaxtyping import Array @@ -361,3 +362,29 @@ def test_jump_at_t1_with_large_t1_in_float32(): saveat=saveat, ) assert sol.ts == jnp.array([t1]) + + +# https://github.com/patrick-kidger/diffrax/issues/713 +def test_t0_at_jump_time(): + jump_time = 0.98 + controller = diffrax.PIDController(rtol=1e-6, atol=1e-6) + controller = diffrax.ClipStepSizeController(controller, jump_ts=[jump_time]) + sol = diffrax.diffeqsolve( + diffrax.ODETerm(lambda t, y, args: jnp.zeros_like(y)), + diffrax.Heun(), + t0=eqxi.prevbefore(jnp.asarray(jump_time)), + t1=1.2, + dt0=None, + y0=jnp.array([0, 0, 0, 0.0]), + stepsize_controller=controller, + event=diffrax.Event( + cond_fn=lambda t, y, args, **kw: jump_time - t, + root_finder=optx.Newton(atol=1e-4, rtol=1e-4), + direction=True, + ), + max_steps=100, + ) + # And in particular not an event. + # What used to happen was something very weird where we'd oscillate across the + # jump time. + assert sol.result == diffrax.RESULTS.successful From fdfecc7a791c874dc627a520d1bedd8394969c72 Mon Sep 17 00:00:00 2001 From: Patrick Kidger <33688385+patrick-kidger@users.noreply.github.com> Date: Fri, 30 Jan 2026 16:31:34 +0100 Subject: [PATCH 21/24] Fix #720; bool event + root find + terminate on first step --- diffrax/_integrate.py | 11 ++++++++--- test/test_event.py | 27 +++++++++++++++++++++++++++ 2 files changed, 35 insertions(+), 3 deletions(-) diff --git a/diffrax/_integrate.py b/diffrax/_integrate.py index 6fc38ce3..3017e30d 100644 --- a/diffrax/_integrate.py +++ b/diffrax/_integrate.py @@ -761,7 +761,11 @@ def _call_real_impl(): _tfinal = _event_root_find.value # TODO: we might need to change the way we evaluate `_yfinal` in order to # get more accurate derivatives? - _yfinal = _interpolator.evaluate(_tfinal) + _yfinal = lax.cond( + final_state.num_steps == 0, + lambda: final_state.y, + lambda: _interpolator.evaluate(_tfinal), + ) _result = RESULTS.where( _event_root_find.result == optx.RESULTS.successful, result, @@ -1323,7 +1327,7 @@ def _allocate_output(subsaveat: SubSaveAt) -> SaveState: event_mask = None else: event_tprev = tprev - event_tnext = tnext + event_tnext = tprev # Fill the dense-info with dummy values on the first step, when we haven't yet # made any steps. # Note that we're threading a needle here! What if we terminate on the very @@ -1334,8 +1338,9 @@ def _allocate_output(subsaveat: SubSaveAt) -> SaveState: # to the end of the interval). # - A floating event can't terminate on the first step (it requires a sign # change). + # c.f. https://github.com/patrick-kidger/diffrax/issues/720 event_dense_info = jtu.tree_map( - lambda x: jnp.empty(x.shape, x.dtype), + lambda x: jnp.zeros(x.shape, x.dtype), dense_info_struct, # pyright: ignore[reportPossiblyUnboundVariable] ) diff --git a/test/test_event.py b/test/test_event.py index 5de0ded6..6a6b97e1 100644 --- a/test/test_event.py +++ b/test/test_event.py @@ -833,3 +833,30 @@ def run(event): t_final, y_final = run(event) assert jnp.allclose(t_final, 10.0) assert jnp.allclose(y_final, jnp.array([10.0, 11.0])) + + +# https://github.com/patrick-kidger/diffrax/issues/720 +def test_boolean_with_root_find_terminating_on_first_step(): + controller = diffrax.PIDController(rtol=1e-6, atol=1e-6) + steady_state_event = diffrax.steady_state_event(rtol=1e-6, atol=1e-6) + root_finder = optx.Newton(atol=1e-4, rtol=1e-4) + + sol = diffrax.diffeqsolve( + diffrax.ODETerm(lambda t, y, args: jnp.zeros_like(y)), + diffrax.Kvaerno5(), + t0=0.0, + t1=1.2, + dt0=None, + y0=jnp.array([10.0]), + stepsize_controller=controller, + event=diffrax.Event( + cond_fn=steady_state_event, + root_finder=root_finder, + ), + saveat=diffrax.SaveAt(t1=True), + max_steps=100, + ) + assert sol.ts is not None + assert sol.ys is not None + assert jnp.allclose(sol.ts, jnp.array([0.0])) + assert jnp.allclose(sol.ys, jnp.array([[10.0]])) From 550c20272360e6bdec1be80b997135357d4dd000 Mon Sep 17 00:00:00 2001 From: andyElking Date: Mon, 29 Jul 2024 17:24:42 +0100 Subject: [PATCH 22/24] Added Advanced SDE example and a table of SRKs --- docs/api/solvers/sde_solvers.md | 5 +- docs/devdocs/SDE_solver_table.md | 46 +++ docs/examples/sde_example.ipynb | 578 +++++++++++++++++++++++++++++++ mkdocs.yml | 2 + 4 files changed, 630 insertions(+), 1 deletion(-) create mode 100644 docs/devdocs/SDE_solver_table.md create mode 100644 docs/examples/sde_example.ipynb diff --git a/docs/api/solvers/sde_solvers.md b/docs/api/solvers/sde_solvers.md index e307c195..9ec5b025 100644 --- a/docs/api/solvers/sde_solvers.md +++ b/docs/api/solvers/sde_solvers.md @@ -1,6 +1,9 @@ # SDE solvers -See also [How to choose a solver](../../usage/how-to-choose-a-solver.md#stochastic-differential-equations). +See also [How to choose a solver](../../usage/how-to-choose-a-solver.md#stochastic-differential-equations) +and [Advanced SDE example](../../examples/sde_example.ipynb) which gives a walkthrough of how to simulate SDEs +and how to perform optimisation with respect to SDE parameters. +For a table of all SDE solvers and their properties see [SDE solver table](../../devdocs/SDE_solver_table.md). !!! info "Term structure" diff --git a/docs/devdocs/SDE_solver_table.md b/docs/devdocs/SDE_solver_table.md new file mode 100644 index 00000000..d1d390eb --- /dev/null +++ b/docs/devdocs/SDE_solver_table.md @@ -0,0 +1,46 @@ +# SDE solver table + +For an explanation of the terms in the table, see [how to choose a solver](../usage/how-to-choose-a-solver.md#stochastic-differential-equations). + +This table is included as a reference that we think *should* be correct, but if you're going to use any of this information in a load-bearing way (for example a publication) then please double-check the information! It's a big table, we might have gotten something wrong. + +``` ++----------------+-------+------------+------------------------------------+-------------------+----------------+------------------------------------------+ +| | SDE | Lévy | Strong/weak order (per noise type) | VF evaluations | Embedded error | Recommended for | +| | type | area +----------+--------------+----------+-------+-----------+ estimation | (and other notes) | +| | | | General | Commutative | Additive | Drift | Diffusion | | | ++----------------+-------+------------+----------+--------------+----------+-------+-----------+----------------+------------------------------------------+ +| Euler | Itô | BM only | 0.5/1.0 | 0.5/1.0 | 1.0/1.0 | 1 | 1 | No | Itô SDEs when a cheap solver is needed. | ++----------------+-------+------------+----------+--------------+----------+-------+-----------+----------------+------------------------------------------+ +| Heun | Strat | BM only | 0.5/1.0 | 1.0/1.0 | 1.0/1.0 | 2 | 2 | Yes | Standard solver for Stratonovich SDEs. | ++----------------+-------+------------+----------+--------------+----------+-------+-----------+----------------+------------------------------------------+ +| EulerHeun | Strat | BM only | 0.5/1.0 | 0.5/1.0 | 1.0/1.0 | 1 | 2 | No | Stratonovich SDEs with expensive drift. | ++----------------+-------+------------+----------+--------------+----------+-------+-----------+----------------+------------------------------------------+ +| ItoMilstein | Itô | BM only | 0.5/1.0 | 1.0/1.0 | 1.0/1.0 | 1 | 1 | No | Better than Euler for Itô SDEs, but | +| | | | | | | | | | comuptes the derivative of diffusion VF. | ++----------------+-------+------------+----------+--------------+----------+-------+-----------+----------------+------------------------------------------+ +| Stratonovich | Strat | BM only | 0.5/1.0 | 1.0/1.0 | 1.0/1.0 | 1 | 1 | No | For commutative Stratonovich SDEs when | +| Milstein | | | | | | | | | space-time Lévy area is not available. | +| | | | | | | | | | Computes derivative of diffusion VF. | ++----------------+-------+------------+----------+--------------+----------+-------+-----------+----------------+------------------------------------------+ +| ReversibleHeun | Strat | BM only | 0.5/1.0 | 1.0/1.0 | 1.0/1.0 | 2 | 2 | Yes | When a reversible solver is needed. | ++----------------+-------+------------+----------+--------------+----------+-------+-----------+----------------+------------------------------------------+ +| Midpoint | Strat | BM only | 0.5/1.0 | 1.0/1.0 | 1.0/1.0 | 2 | 2 | Yes | Usually Heun should be preferred. | ++----------------+-------+------------+----------+--------------+----------+-------+-----------+----------------+------------------------------------------+ +| Ralston | Strat | BM only | 0.5/1.0 | 1.0/1.0 | 1.0/1.0 | 2 | 2 | Yes | Usually Heun should be preferred. | ++----------------+-------+------------+----------+--------------+----------+-------+-----------+----------------+------------------------------------------+ +| ShARK | Strat | space-time | / | / | 1.5/2.0 | 2 | 2 | Yes | Additive noise SDEs. | ++----------------+-------+------------+----------+--------------+----------+-------+-----------+----------------+------------------------------------------+ +| SRA1 | Strat | space-time | / | / | 1.5/2.0 | 2 | 2 | Yes | Only slightly worse than ShARK. | ++----------------+-------+------------+----------+--------------+----------+-------+-----------+----------------+------------------------------------------+ +| SEA | Strat | space-time | / | / | 1.0/1.0 | 1 | 1 | No | Cheap solver for additive noise SDEs. | ++----------------+-------+------------+----------+--------------+----------+-------+-----------+----------------+------------------------------------------+ +| SPaRK | Strat | space-time | 0.5/1.0 | 1.0/1.0 | 1.5/2.0 | 3 | 3 | Yes | General SDEs when embedded error | +| | | | | | | | | | estimation is needed. | ++----------------+-------+------------+----------+--------------+----------+-------+-----------+----------------+------------------------------------------+ +| GeneralShARK | Strat | space-time | 0.5/1.0 | 1.0/1.0 | 1.5/2.0 | 2 | 3 | No | General SDEs when embedded error | +| | | | | | | | | | estimaiton is not needed. | ++----------------+-------+------------+----------+--------------+----------+-------+-----------+----------------+------------------------------------------+ +| SlowRK | Strat | space-time | 0.5/1.0 | 1.5/2.0 | 1.5/2.0 | 2 | 5 | No | Commutative noise SDEs. | ++----------------+-------+------------+----------+--------------+----------+-------+-----------+----------------+------------------------------------------+ +``` \ No newline at end of file diff --git a/docs/examples/sde_example.ipynb b/docs/examples/sde_example.ipynb new file mode 100644 index 00000000..c00e56b2 --- /dev/null +++ b/docs/examples/sde_example.ipynb @@ -0,0 +1,578 @@ +{ + "cells": [ + { + "cell_type": "code", + "execution_count": 1, + "id": "initial_id", + "metadata": { + "ExecuteTime": { + "end_time": "2025-10-09T01:03:28.847571Z", + "start_time": "2025-10-09T01:03:28.174400Z" + }, + "collapsed": true + }, + "outputs": [ + { + "name": "stdout", + "output_type": "stream", + "text": [ + "env: JAX_PLATFORM_NAME=cuda\n" + ] + } + ], + "source": [ + "%env JAX_PLATFORM_NAME=cuda\n", + "\n", + "from warnings import simplefilter\n", + "\n", + "\n", + "simplefilter(\"ignore\", category=FutureWarning)\n", + "\n", + "from functools import partial\n", + "\n", + "import diffrax\n", + "import equinox as eqx\n", + "import jax\n", + "import jax.numpy as jnp\n", + "import jax.random as jr\n", + "import matplotlib.pyplot as plt\n", + "import optax\n", + "from jaxtyping import Array\n", + "\n", + "\n", + "jax.config.update(\"jax_enable_x64\", True)\n", + "jnp.set_printoptions(precision=4, suppress=True)" + ] + }, + { + "cell_type": "markdown", + "id": "86d4e8b062a81d7e", + "metadata": {}, + "source": [ + "# Advanced SDE example\n", + "\n", + "We will be simulating a Stratonovich SDE of the form:\n", + "\n", + "$$\n", + " dy(t) = f(y(t), t) dt + g(y(t), t) \\circ dw(t), \n", + "$$\n", + "\n", + "where $t \\in [0, T]$, $y(t) \\in \\mathbb{R}^e$, and $w$ is a standard Brownian motion on $\\mathbb{R}^d$. We refer to $f: \\mathbb{R}^e \\times [0, T] \\to \\mathbb{R}^e$ as the drift vector field and $g: \\mathbb{R}^e \\times [0, T] \\to \\mathbb{R}^{e \\times d}$ is the diffusion matrix field. The Stratonovich integral is denoted by $\\circ$.\n", + "\n", + "Our SDE will have the following drift and diffusion terms:\n", + "\n", + "\\begin{align*}\n", + " f(y(t), t) &= \\alpha - \\beta y(t), \\\\\n", + " g(y(t), t) &= \\gamma \\begin{bmatrix} \\Vert y(t) \\Vert_2 & 0 \\\\ 0 & y_1(t) \\\\ 0 & 10t \\end{bmatrix},\n", + "\\end{align*}\n", + "\n", + "where $\\alpha, \\gamma \\in \\mathbb{R}^3$ and $\\beta \\in \\mathbb{R}_{\\geq 0}$ are some parameters.\n", + "\n", + "Let's write the SDE in the form that Diffrax expects:" + ] + }, + { + "cell_type": "code", + "execution_count": 2, + "id": "ba23e9cc0370fbac", + "metadata": { + "ExecuteTime": { + "end_time": "2025-10-09T01:03:29.331509Z", + "start_time": "2025-10-09T01:03:28.853859Z" + } + }, + "outputs": [], + "source": [ + "# Drift VF (e = 3)\n", + "def f(t, y, args):\n", + " alpha, beta, gamma = args\n", + " beta = jnp.abs(beta)\n", + " assert alpha.shape == (3,)\n", + " return jnp.array(alpha - beta * y, dtype=y.dtype)\n", + "\n", + "\n", + "# Diffusion matrix field (e = 3, d = 2)\n", + "def g(t, y, args):\n", + " alpha, beta, gamma = args\n", + " assert gamma.shape == y.shape == (3,)\n", + " gamma = jnp.reshape(gamma, (3, 1))\n", + " out = gamma * jnp.array(\n", + " [[jnp.sqrt(jnp.sum(y**2)), 0.0], [0.0, 3 * y[0]], [0.0, 20 * t]], dtype=y.dtype\n", + " )\n", + " return out\n", + "\n", + "\n", + "# Initial condition\n", + "y0 = jnp.array([1.0, 1.0, 1.0])\n", + "\n", + "# Args\n", + "alpha = 0.5 * jnp.ones((3,))\n", + "beta = 1.0\n", + "gamma = jnp.ones((3,))\n", + "args = (alpha, beta, gamma)\n", + "\n", + "# Time domain\n", + "t0 = 0.0\n", + "t1 = 2.0\n", + "dt0 = 2**-9" + ] + }, + { + "cell_type": "markdown", + "id": "ef2ff90865907b7d", + "metadata": {}, + "source": [ + "## Brownian motion and its Levy area\n", + "\n", + "Different solvers require different information about the Brownian motion. For example, the `SPaRK` solver requires access to the space-time Levy area of the Brownian motion. The required Levy area for each solver is documented in the table at the end of this notebook, or can be checked via `solver.minimal_levy_area`.\n", + " \n", + "We will use the `VirtualBrownianTree` class to generate the Brownian motion and its Levy area." + ] + }, + { + "cell_type": "code", + "execution_count": 3, + "id": "4110735158215acc", + "metadata": { + "ExecuteTime": { + "end_time": "2025-10-09T01:03:29.483519Z", + "start_time": "2025-10-09T01:03:29.337297Z" + } + }, + "outputs": [ + { + "name": "stdout", + "output_type": "stream", + "text": [ + "Minimal levy area for SPaRK: .\n" + ] + } + ], + "source": [ + "# check minimal levy area\n", + "solver = diffrax.SPaRK()\n", + "print(f\"Minimal levy area for SPaRK: {solver.minimal_levy_area}.\")\n", + "\n", + "# Brownian motion\n", + "key = jr.key(0)\n", + "bm_tol = 2**-13\n", + "bm_shape = (2,)\n", + "bm = diffrax.VirtualBrownianTree(\n", + " t0, t1, bm_tol, bm_shape, key, levy_area=diffrax.SpaceTimeLevyArea\n", + ")\n", + "\n", + "# Defining the terms of the SDE\n", + "ode_term = diffrax.ODETerm(f)\n", + "diffusion_term = diffrax.ControlTerm(g, bm) # Note that the BM is baked into the term\n", + "terms = diffrax.MultiTerm(ode_term, diffusion_term)" + ] + }, + { + "cell_type": "markdown", + "id": "e71db03c5257bd46", + "metadata": {}, + "source": [ + "### Using `diffrax.diffeqsolve` to solve the SDE\n", + "\n", + "We will first use constant steps of size $h = 2^{-9}$ to solve the SDE. It is very important to have $h > \\mathtt{bm\\_tol}$, where $\\mathtt{bm\\_tol}$ is the tolerance of the Brownian motion. This is important because the output distribution of the VirtualBrownianTree is precise as long as the times that we sample it at are at least $\\mathtt{bm\\_tol}$ apart. For more details see the [Single-seed Brownian Motion paper](https://arxiv.org/abs/2405.06464).\n", + "\n", + " We will use the SPaRK solver to solve the SDE. SPaRK is a stochastic Runge-Kutta method that requires access to space-time Levy area." + ] + }, + { + "cell_type": "code", + "execution_count": 4, + "id": "8a969e1b9bd9f09", + "metadata": { + "ExecuteTime": { + "end_time": "2025-10-09T01:03:35.618860Z", + "start_time": "2025-10-09T01:03:29.493121Z" + } + }, + "outputs": [ + { + "data": { + "image/png": "iVBORw0KGgoAAAANSUhEUgAAAhAAAAKqCAYAAABrUWeaAAAAOXRFWHRTb2Z0d2FyZQBNYXRwbG90bGliIHZlcnNpb24zLjkuMiwgaHR0cHM6Ly9tYXRwbG90bGliLm9yZy8hTgPZAAAACXBIWXMAAA9hAAAPYQGoP6dpAAEAAElEQVR4nOydd3wT5R/HP9npTHdLF2WXVfYWmTJFQMQFCAi4wIUTJ07coP5QFAciggqKgmyRIZRZhswyS0v33k3S5H5/XO6SSy5p0pWO7/v1yqt3zz13eVJK7nPfKWEYhgFBEARBEIQLSN29AIIgCIIgGh8kIAiCIAiCcBkSEARBEARBuAwJCIIgCIIgXIYEBEEQBEEQLkMCgiAIgiAIlyEBQRAEQRCEy5CAIAiCIAjCZUhAEARBEAThMiQgCIJoMCxevBgSiaTWrxsTE4NZs2bV+nUJojlDAoIgGglnzpzBXXfdhZYtW0KtViMiIgK33XYbPv/8c8G8mJgYSCQSSCQSSKVS+Pn5oWvXrnjooYdw5MgR0Wtz88VejzzySH18vBoTHx+PxYsXo6CgwN1LIYhmgdzdCyAIomri4+MxbNgwREdHY968eQgLC0NKSgoOHz6MTz/9FI8//rhgfvfu3fHMM88AAIqLi3HhwgWsX78eK1euxNNPP41PPvnE5j1uu+02PPDAAzbj7du3r5sPVcvEx8fjjTfewKxZs+Dn5yc4lpiYCKmUnpcIojYhAUEQjYB33nkHGo0Gx44ds7k5ZmVl2cyPiIjA9OnTBWPvv/8+7r//fixduhTt2rXDo48+Kjjevn17m3OaCiqVyt1LIIgmB0lygmgEXL16FZ07d7YRDwAQEhLi1DU8PDzw448/IiAgAO+88w5qqxFvcXExnnrqKcTExEClUiEkJAS33XYbTpw4IZi3fv169OrVCx4eHggKCsL06dORmprq8NpJSUmQSCRYtWqVzTGJRILFixcDYGMnnnvuOQBAq1atePdLUlISAPEYiGvXrmHq1KkICAiAp6cn+vfvjy1btgjm7N27FxKJBL/++iveeecdREZGQq1WY8SIEbhy5YrzvySCaIKQBYIgGgEtW7bEoUOHcPbsWXTp0qXa1/H29sbkyZPx7bff4vz58+jcuTN/rKKiAjk5OTbn+Pr6QqlU2r3mI488gg0bNmDBggXo1KkTcnNzceDAAVy4cAE9e/YEAKxatQqzZ89Gnz59sGTJEmRmZuLTTz/FwYMHcfLkSVFh5Ap33nknLl26hHXr1mHp0qUICgoCAAQHB4vOz8zMxMCBA1FWVoYnnngCgYGB+OGHH3DHHXdgw4YNmDx5smD+e++9B6lUimeffRaFhYX44IMPMG3aNLsxJQTRLGAIgmjw7Ny5k5HJZIxMJmMGDBjAPP/888yOHTsYnU5nM7dly5bM+PHj7V5r6dKlDADmzz//5McA2H2tW7fO4do0Gg0zf/58u8d1Oh0TEhLCdOnShSkvL+fH//rrLwYA89prr/Fjr7/+OmP5tXT9+nUGAPP999/bXBcA8/rrr/P7H374IQOAuX79us3cli1bMjNnzuT3n3rqKQYA8++///JjxcXFTKtWrZiYmBjGYDAwDMMwe/bsYQAwHTt2ZLRaLT/3008/ZQAwZ86csfu5CaKpQy4MgmgE3HbbbTh06BDuuOMOnD59Gh988AFGjx6NiIgIbNq0yaVreXt7A2BdD5ZMnDgRu3btsnkNGzbM4fX8/Pxw5MgRpKWliR4/fvw4srKy8Nhjj0GtVvPj48ePR2xsrI3boD7YunUr+vbti1tuuYUf8/b2xkMPPYSkpCScP39eMH/27NkCK8zgwYMBsG4QgmiukAuDIBoJffr0we+//w6dTofTp09j48aNWLp0Ke666y6cOnUKnTp1cuo6JSUlAAAfHx/BeGRkJEaOHOnyuj744APMnDkTUVFR6NWrF8aNG4cHHngArVu3BgDcuHEDANChQwebc2NjY3HgwAGX37Om3LhxA/369bMZ79ixI3/c0lUUHR0tmOfv7w8AyM/Pr8NVEkTDhiwQBNHIUCqV6NOnD9599118+eWX0Ov1WL9+vdPnnz17FgDQtm3bWlnP3XffjWvXruHzzz9HeHg4PvzwQ3Tu3Bnbtm2r8bXtFZUyGAw1vrYryGQy0XGmlgJRCaIxQgKCIBoxvXv3BgCkp6c7Nb+kpAQbN25EVFQU/7RdG7Ro0QKPPfYY/vjjD1y/fh2BgYF45513ALABoABbi8GaxMRE/rgY3JO+dXEozqphiSsVLFu2bCm6nosXLwrWTBCEfUhAEEQjYM+ePaJPu1u3bgUg7h6wpry8HDNmzEBeXh5efvnlWikZbTAYUFhYKBgLCQlBeHg4tFotAFbkhISEYMWKFfwYAGzbtg0XLlzA+PHj7V7f19cXQUFB2L9/v2D8iy++sJnr5eUFwFZsiDFu3DgcPXoUhw4d4sdKS0vx9ddfIyYmxml3EEE0ZygGgiAaAY8//jjKysowefJkxMbGQqfTIT4+Hr/88gtiYmIwe/ZswfzU1FSsWbMGAGt1OH/+PNavX4+MjAw888wzePjhh23e49KlS/w5loSGhuK2224TXVdxcTEiIyNx1113oVu3bvD29sbff/+NY8eO4eOPPwYAKBQKvP/++5g9ezaGDBmC++67j0/jjImJwdNPP+3ws8+dOxfvvfce5s6di969e2P//v24dOmSzbxevXoBAF5++WXce++9UCgUmDBhAi8sLHnxxRexbt06jB07Fk888QQCAgLwww8/4Pr16/jtt9+oaiVBOIO700AIgqiabdu2MQ8++CATGxvLeHt7M0qlkmnbti3z+OOPM5mZmYK5LVu25FMwJRIJ4+vry3Tu3JmZN28ec+TIEdHrw0Ea55AhQ+yuS6vVMs899xzTrVs3xsfHh/Hy8mK6devGfPHFFzZzf/nlF6ZHjx6MSqViAgICmGnTpjE3b94UzLFO42QYhikrK2PmzJnDaDQaxsfHh7n77ruZrKwsmzROhmGYt956i4mIiGCkUqkgpdM6jZNhGObq1avMXXfdxfj5+TFqtZrp27cv89dffwnmcGmc69evF4w7Si8liOaChGEoCoggCIIgCNcgOx1BEARBEC5DAoIgCIIgCJchAUEQBEEQhMuQgCAIgiAIwmVIQBAEQRAE4TIkIAiCIAiCcJkmV0jKaDQiLS0NPj4+tVJpjyAIgiCaCwzDoLi4GOHh4VUWVGtyAiItLQ1RUVHuXgZBEARBNFpSUlIQGRnpcE6TExBci+KUlBT4+vq6eTUEQRAE0XgoKipCVFQUfy91RJMTEJzbwtfXlwQEQRAEQVQDZ0IAKIiSIAiCIAiXIQFBEARBEITLkIAgCIIgCMJlmlwMhLMYDAbo9Xp3L6NBolQqq0zfIQiCIJo3zU5AMAyDjIwMFBQUuHspDRapVIpWrVpBqVS6eykEQRBEA6XZCQhOPISEhMDT05OKTVnBFeJKT09HdHQ0/X4IgiAIUZqVgDAYDLx4CAwMdPdyGizBwcFIS0tDZWUlFAqFu5dDEARBNECalaObi3nw9PR080oaNpzrwmAwuHklBEEQREOlWQkIDjLLO4Z+PwRBEA0LhmHcvQQbmqWAIAiCIIjGwh9X/sCgdYNwLOOYu5cigAQEQRAEQTRgXj34Kor1xXh+//PuXooAEhBNnPT0dNx///1o3749pFIpnnrqKXcviSAIgqgGOeU57l6CABIQTRytVovg4GC88sor6Natm7uXQxAEQdSA4xnH3b0EHhIQjYDVq1cjMDAQWq1WMD5p0iTMmDHD4bkxMTH49NNP8cADD0Cj0dTlMgmCIIhapkhXJNjfen2rm1ZiS7OqAyEGwzAo17snXdFDIXMq42Hq1Kl44oknsGnTJkydOhUAkJWVhS1btmDnzp11vUyCIAjCTWy+ulmwn1+R76aV2NLsBUS53oBOr+1wy3uff3M0PJVV/xN4eHjg/vvvx/fff88LiDVr1iA6OhpDhw6t41USBEEQ7iK3PFewX6AtcM9CRCAXRiNh3rx52LlzJ1JTUwEAq1atwqxZs6hmA0EQRBOmUFsIAIgLigMAlOpL3bkcAc3eAuGhkOH8m6Pd9t7O0qNHD3Tr1g2rV6/GqFGjcO7cOWzZsqUOV0cQBEG4m0IdKyBaeLfAfzn/oURf4uYVmWn2AkIikTjlRmgIzJ07F8uWLUNqaipGjhyJqKgody+JIAiCqEM4C0S4dziAhmWBIBdGI+L+++/HzZs3sXLlSjz44INOn3fq1CmcOnUKJSUlyM7OxqlTp3D+/Pk6XClBEARRG3BZGOFeDU9ANI5HbwIAoNFoMGXKFGzZsgWTJk1y+rwePXrw2wkJCVi7di1atmyJpKSk2l8kQRAEUWtYWyC0Bi10Bh2UMqU7lwWALBCNjtTUVEybNg0qlcrpcxiGsXmReCAIgmj4FGmFFggA+Pbst+5ajgASEI2E/Px8bNy4EXv37sX8+fPdvRyCIAiijsgqy8L2pO0o1ZeiWF8MAAjwCOCPb7nWMALoyYXRSOjRowfy8/Px/vvvo0OHDvx4586dcePGDdFzvvrqK0ybNq2+lkgQBEHUEIZhMGL9CJtxH6UP3h70Nl45+AqKdcVuWJktJCAaCfZcDlu3boVerxc9FhoaWocrIgiCIGqbpKIk0XGFVIF+LfoBsC1v7S5IQDRyWrZs6e4lEARBELWEdeVJS9QyNQCg0liJSmMl5FL33sIpBoIgCIIgGghc4SgxVHJz8LzWoLU7r74gAUEQBEEQDQQubdOSdePXAQBUMrOAqKisqLc12cOtAuK9996DRCLBU089xY9VVFRg/vz5CAwMhLe3N6ZMmYLMzEz3LZIgCIJwH5VaQNswggbrA7FmWbEBsQAAqUTKi4hmbYE4duwYvvrqK8TFxQnGn376aWzevBnr16/Hvn37kJaWhjvvvNNNqyQIgiDcyk9TgWVdgZJsd6+kXhCzQFjGOnACosLQTC0QJSUlmDZtGlauXAl/f39+vLCwEN9++y0++eQTDB8+HL169cL333+P+Ph4HD582B1LJQiCINxFeT5wfR/787+fgX/eBnIuu3tVdUpKcYrD41wgZbN1YcyfPx/jx4/HyJEjBeMJCQnQ6/WC8djYWERHR+PQoUP1vUyCIAjCnaSfNm/vfAXY/yGwfrb71lOHcFWCz+c67lOklrMCQmvQ4odzP+DTE58iuSi5PpZoQ73ngPz88884ceIEjh07ZnMsIyMDSqUSfn5+gvHQ0FBkZGSIXk+r1UKrNfuCiooaRn5sQ+H333/Hl19+iVOnTkGr1aJz585YvHgxRo92TwtzgiAIpynJsh3LPAMwDCCR1P966giGYfDgjgchkUiQWpLqcC6XiVFRWYHfL/+Oa4XXMDB8IKJ9o+tjqQLq1QKRkpKCJ598Ej/99BPUanWtXHPJkiXQaDT8i1pcC9m/fz9uu+02bN26FQkJCRg2bBgmTJiAkydPuntpBEEQjtFaPBCq/czbxem1c/3yfCDN/d+FmWWZOJ55HMcybB+sQzxDBPuWLgy9kS0iqJAq6n6RItSrgEhISEBWVhZ69uwJuVwOuVyOffv24bPPPoNcLkdoaCh0Oh0KCgoE52VmZiIsLEz0mosWLUJhYSH/Sklx7D9qjKxevRqBgYECSwsATJo0CTNmzHB47rJly/D888+jT58+aNeuHd599120a9cOmzdvrsslEwRB1Bwu+6L7NODFG4DG9JRdUM3v+ZIsYOerwM3j7P7mJ4GvhwJnNtR4qTVBZ9DZPbZy1ErBvqULgxcQsmYgIEaMGIEzZ87g1KlT/Kt3796YNm0av61QKLB7927+nMTERCQnJ2PAgAGi11SpVPD19RW8XIJhAF2pe14M49QSp06dCoPBgE2bNvFjWVlZ2LJlCx588EGXPq7RaERxcTECAgKqnkwQBOFOOAGh8mF/+pkszIXVFBAJq4D4z4BvRgBlecD5P9nxrc/VaJk1pVRfavdYa01rwb5lFgYnPNxlgajXGAgfHx906dJFMObl5YXAwEB+fM6cOVi4cCECAgLg6+uLxx9/HAMGDED//v3rZlH6MuDd8Krn1QUvpQFKryqneXh44P7778f333+PqVOnAgDWrFmD6OhoDB061KW3/Oijj1BSUoK77767OismCIKoP6wFhKaGAiLvunn73Ebzdnle9a5XQxIyE/DWobcwMGKg0+dwLgxtpdkCoZQq62R9VdHgemEsXboUUqkUU6ZMgVarxejRo/HFF1+4e1luZ968eejTpw9SU1MRERGBVatWYdasWZC4EEi0du1avPHGG/jzzz8REhJS9QkEQRDuRFvC/lR6sz+9gtif5fmuX6vwJnB6rXn/+n7h8YpCQK1x/brVpFhXjFnbZwEArhZedfo8PojSUAG9wb0uDLcLiL179wr21Wo1li9fjuXLl9fPAhSerCXAHSg8nZ7ao0cPdOvWDatXr8aoUaNw7tw5bNnifE/4n3/+GXPnzsX69ett0mcJgiAaJFwQJWeBUJlc1BX2+0XYJXGbcD/liHA/7xoQ3sP161aTmdtnVuu8hhRE6XYB4XYkEqfcCA2BuXPnYtmyZUhNTcXIkSOdzjhZt24dHnzwQfz8888YP358Ha+SIAiiluAFhEk4cBaCimqk61sHKlpncvzxGPDQXsCiYVVdcjnffkGsezrcg9iAWPQJ62NzjAuiLKssg4ExAHCfC4OaaTUi7r//fty8eRMrV650Onhy7dq1eOCBB/Dxxx+jX79+yMjIQEZGBgoLq6HgCYIg6pMyU2yCpynomxcQLn5/5V4FdrzEbveYLj4n6zzwaTeng9trAlPFe/gqfXFX+7vQ0relzTEuiDI+LZ4faxZZGETN0Gg0mDJlCry9vTFp0iSnzvn6669RWVmJ+fPno0WLFvzrySefrNvFEgRB1JTSHPanZyD7U22yRFzdDWx/CTDFAFTJWougcU202aJhTXE6cHlX9dbqAsV6x83BPB24tzkLxMW8i/wYWSAIp0hNTcW0adOgUjlnZtu7dy9fItXytWrVqrpdKEEQzYOyPMBoqP3rMgxQlstuc8GTlkGOh5cDf86v+jqXdwG5V8z7XoGAd6h5/5aFwvmW5bPriMxSxx2mPeUOBITMtgijZbOt+oQERCMhPz8fGzduxN69ezF/vhP/aQiCIOqanMvAh22B3+fV/rW1RYApSJC3QPi0EM757xcgO5Hdzr0K/PU0m23BkXUR+Oku4TnhPQAfi8KELeKA+34x7+ck1s76HWBpPRDDkQWCc2FY4ko2Xm1CAqKR0KNHD8yaNQvvv/8+OnTowI937twZ3t7eoq+ffvrJjSsmCKLJc+QrgDEAZ39z/dz8G8Dmp4D0/8SPc9YHhReg8GC3A9sA7UYJ5x3/jv256Ql2e80U87ETq83bd64EnjoLRPQSWiB8I4EOY4Bu97P7OZeBojTn3SPVoKqGWTKJzO4xD7lHbS+n2lAWRiMhKSlJdHzr1q3Q68X/0ENDQ0XHCYIgagWZhe9dX26+0VvDMEDyYdYFkXWetQJ83pM9dvY3YJFIYahSk4DgrA8c964DTq5m3SZbnwVOrgEGPg7cOMAez77IWiMC2wDZF9ixO/4HxFnEQViuUxPJ/rzlKbZORPop4JOOwIAFwOh3nPktuExSUZLNWLRPNJKL2a6arlog3AUJiEZOy5a2UboEQRD1gr7MvL3lGdZqMOUbc90Gjn8/Av5527yvtDiuLQKKM4RuBQAoMwVQelkJCJkc6P0gYDQCh79g6zcs7Sycc+VvVkDk32D3/WOEx7nCVID5ff1bsYKIS/c89L86ExA3i2/ajIV4hmDxwMU4lnEMQyKH2D3XS9Fwyg6QC4MgCIKoHiUWwYCnfgIubQeWRAJFphoLhkrg0g6heAAAnVUWAueGsCT+f+xPawsEh1QK3L5MONZhHPvz0g5gsQbIM1V49Ld60Bq8EIi9HZi1xdwWXK4EAtoI5yX8wP688jdQkCy+DhfRG/W4WWIrIF7t/yr6hPXBY90fcxgU2SGgg91j9Q0JCIIgCKJ6WFogLPkklnVbnN0gTKG0RKYC+j7MbudcEh6r1JpdEpYZFNa0HgJM+RYIaA2MeB3o9wg7ftXckBHhPdg4B0u8Q4B7fwJibhGO950rtE5sfoIVImumsF07a4FzOedQaayEr9KcSvr96O/R2q+1g7PMhHmJd6Z2ByQgCIIgiOpRqbV/7Oo/bDCiPaL7A22Gsdu5Vr0gLJteeQY5XkPXu4AnTrJWhVa3ApF9hcfv+Ym1VjhDn7nAiynAqzls8KYlXFBnDeECKHuG9MTQyKGI9olG1+CuLl3j21HfIsI7AgDwXG/3dRKlGAiCIAiielRW2D+WfAhwVJ8gvAcbdwCYYxU4ci3KPN/xufPrkUiAbvcCN4+y+wMfBzQRzp8PmMSGFJj6va31xGgApPYzJJyBa93tr/bHGwPfgJExQubiNfu26IvtU7bXaB21AVkgCIIgiOrhyAKRdYHNzLCHZyDgG85uawvNZasBNoYBAHrOBEI7ubamYIsYAWvXhSu0Hw3c9pZwzJFFxUnKK9nfiafCExKJxGXx0JAgAUEQBEFUD2sLRJsRbEwCwAYdOhIYgW3NpakBYJWp0d+xb4CTP7Lb7ce4vqaQToCp3DM6jHX9fEv6zAUC25n3a6FgVlklGzfSkOo5VBcSEE2cAwcOYNCgQQgMDISHhwdiY2OxdOlSdy+LIIimgLVAGLMECDLdcEuybAXG9N+Bu74HBj1lKw6yzgM5V9h0UA7r7Aln8AwAHtwBzD9avfMtUXoCjx9n+2cArFumUuf4nCrgLBBNQUBQDEQTx8vLCwsWLEBcXBy8vLxw4MABPPzww/Dy8sJDDz3k7uURBNGYsRYIvhGAroTdLs0WZmk8uBOI7sdud7nTPN7lLjZbAwD+10t4Pe9qFsML71698+zxeALwdjC7nXsZCO3seL4DyvVNR0CQBaIRsHr1agQGBkKrFar9SZMmYcaMGQ7P7dGjB+677z507twZMTExmD59OkaPHo1///23LpdMEERzgLNA3PE5W1NB5W3KmpCwJa4LU9njYz8wiwdrJiyzf32PgNpcbfWRK83ZHTXs1snHQDhomNVYaPYCgmEYlOnL3PKqqic8x9SpU2EwGLBp0yZ+LCsrC1u2bMGDDz7o0uc9efIk4uPjMWSI/UpnBEEQTsFZINreZq6pIJOzbgQAKDSVqJY7KL+s8mHjFsRwNv2yPgjrwv489i1b46KaNKUYiGbvwiivLEe/tXaUcR1z5P4jDmuec3h4eOD+++/H999/j6lTpwIA1qxZg+joaAwdOtSp94qMjER2djYqKyuxePFizJ07tyZLJwiiuVGpA1bfwcYBTP8NaD0MMFayx6wFgsqHrZvA1U6o6mYp1kPj1udrvuba5Nbn2IqZhclsgGg14yuaUgxEA5J3hCPmzZuHnTt3IjWVNQmuWrUKs2bNcrqN67///ovjx49jxYoVWLZsGdatW1eXyyUIoqlxdTcrHgBg4yPCAEou64GDq+bIxUM4skAAgNiD1KAnq7fOusI3nK1dAQBpJ6p9Gcs0zsZOs7dAeMg9cOT+I257b2fp0aMHunXrhtWrV2PUqFE4d+4ctmzZ4vT5rVqxBVu6du2KzMxMLF68GPfdd5/LayYIoplSaNG/QSITBlBaCwSlVRVHa4FhjWVXTwCY+AUbT9HQCOoApJ20rZzpAk3JAtHsBYREImk0SnDu3LlYtmwZUlNTMXLkSERFRVXrOkaj0SYgkyAIwiH5SeZt72CzgJDKbaszWn+nKqoQEAaL1MhXcwCZotrLrFMC27I/D/0PuHmMFTrW3UKroCkJCHJhNCLuv/9+3Lx5EytXrnQ6eHL58uXYvHkzLl++jMuXL+Pbb7/FRx99hOnTp9fxagmCaBKc2QB8Nwa4edw8VlEIlBew29bWBrGxqtIxKwrN2w1VPADmKpfl+Wzn0b9fd/kSZXoKoiTcgEajwZQpU7BlyxZMmjTJqXOMRiMWLVqE69evQy6Xo02bNnj//ffx8MMP1+1iCYJoGvw2x3asohD4coBpRyQOy7KjpdoPCOno+D1Kc6q7uvolur9w/+SPwC1PA4FtxOdbYWSMqDCwlpvGYvl2BAmIRkZqaiqmTZsGlaqKoCQTjz/+OB5//PE6XhVBEM0KS4tBRYHtcUsLRG8nrKVljURAeIcAnSYC5/80j21fBEz71anTKyziRpqCBYJcGI2E/Px8bNy4EXv37sX8+fPdvRyCIAj7FGeYt52p2th9Gvuz3ai6WU9tcvdq4IUbwO3L2P3LO4ALfzl1KlcDQgIJ1LIq4kIaAfUqIL788kvExcXB19cXvr6+GDBgALZt28Yfr6iowPz58xEYGAhvb29MmTIFmZmZ9bnEBkuPHj0wa9YsvP/+++jQwdxtrnPnzvD29hZ9/fTTT25cMUEQjRqjEbi2t3rndpzA/mw/Fuh8p+O5ADD6HbYJF9eIq6Hj4Qf0ng30Nz3M/TINiP8cyDjj8DTLAEpnU/AbMvXqwoiMjMR7772Hdu3agWEY/PDDD5g4cSJOnjyJzp074+mnn8aWLVuwfv16aDQaLFiwAHfeeScOHjxYn8tskCQlJYmOb926FXq9XvRYaGg168gTBEGc+AH46ynb8ZjBQFIVpfDj7gEierGNtZy5USq9gK53VWuZbmXAY2z3UIMW2PkK4NMCWHjB7mc+kHoAgNkS0dipVwExYcIEwf4777yDL7/8EocPH0ZkZCS+/fZbrF27FsOHDwcAfP/99+jYsSMOHz6M/v37i12y2dOyZQ27zREEQYhxzI41oOfMqgWEVAoEt6/9NTU0NJHAnV8Dh5YDN48CxelAzmW7n/3dI+/W8wLrFrfFQBgMBvz8888oLS3FgAEDkJCQAL1ej5EjR/JzYmNjER0djUOHDrlrmQRBEM0TbZFwf+yHwOSvbTMRuPiF5krnScDcXUC0KSvFTpXKpMIkfruNxrmsjYZOvWdhnDlzBgMGDEBFRQW8vb2xceNGdOrUCadOnYJSqYSfn59gfmhoKDIyMsQvBkCr1QqKIhUVFdmdy+FsE6vmCv1+CKKZo68ACm6Y9yUyoO881jSvKzWPB3UAxn9c/+triATHsqW+cy4BSQeADXOAIc8Bfdi+Q79f/h0AEKAOwI/jfnTnSmuNerdAdOjQAadOncKRI0fw6KOPYubMmTh//ny1r7dkyRJoNBr+5ag6o0LBFigpK2sa/qe6Qqdjq8LJZLIqZhIE0SSxDgb0CjL79S3rF7S6VbwRVnOEKzL178fAqvFASQaw5Rkg6wJw/Huc/281AOCJHk/AR+njxoXWHvVugVAqlWjbli0H2qtXLxw7dgyffvop7rnnHuh0OhQUFAisEJmZmQgLC7N7vUWLFmHhwoX8flFRkV0RIZPJ4Ofnh6ysLACAp6dnk4iErU2MRiOys7Ph6ekJuZzKhBBEs0Rv9ZAVHGvetvzOtCxB3dxpPxrY/RagLxWOn/4ZFfGf4mx0BAApYgNjRU9vjLj9DsH1ZejVqxcUCgV2796NKVOmAAASExORnJyMAQMG2D1fpVI5XVQJAC9GOBFB2CKVShEdHU3iiiCaK5aNsgD7lSQN4hlgzZKA1sC8f4BvbwMgAVrfClzYDKQm4LhahVKpFGFQoGNAFVU5GxH1KiAWLVqEsWPHIjo6GsXFxVi7di327t2LHTt2QKPRYM6cOVi4cCECAgLg6+uLxx9/HAMGDKjVDAyJRIIWLVogJCTEbvpjc0epVEIqpRpjBNFs0ZcL9/1jhPv9HgUSVgGDn6mvFTUOQmKBRSnsdsYZ4MJmHM49h0fDQgAAMRIVpJKm891arwIiKysLDzzwANLT06HRaBAXF4cdO3bgtttuAwAsXboUUqkUU6ZMgVarxejRo/HFF1/UyVpkMhn5+AmCIMSoSkCMfQ8YubjqLpvNmaD2gESKT33MzcHUTUg8APUsIL791nGVMbVajeXLl2P58uX1tCKCIAjChkoLAREWB7QZbjuHxINj5CowAa1xVmWOE1FLmtZDa9OSQwRBEETN0ZtiILrcBTy8H5A7H2dGmCnpeLtgXybWubQRQwKCIAiCEMJZIBQezpWirmOMjFGwn1eRhx/O/YDc8lw3rcg58sK7CvaNDeB3WZuQgCAIgiCE6C0EhJtZHL8YvX7shb+umTtePr//eXx0/CMs3LvQwZnuJ0+mFOwb7cxrrJCAIAiCIIRwAkLu3jgHg9GA3y7/hkqmEjuSdvDjR9KPAABOZImXjW4oJBuFNSFIQBAEQRBNG64OhGXVSTeQUWa/jUFjYGvKXsG+sWl5MNxfSIogCIJoIBj0gERq4cJwrwXiRpG5H0e5aU0NPe7BkiuFVwT7Ta3NEAkIgiAIAqjUAcv7At4hgIc/O6Zyb8+GgooCfrtEXwIAGPrrUPcsxkUqjZXIKc8RjBkZg5tWUzeQgCAIgiDYyon519kXh0rjvvUAKLeoR1GiLxFYJBo6OeU5MDJGyCFFpSn6oa2kaaXDUgwEQRAEARhFSvu72QIhEBC6EuRV5NX7GoyMEd+d/Q6nsk45fc6htEN458g7AIBgr1D81Oo+zC4owjwE1NEq3QNZIAiCIAhAV2I7pvat/3VYYG2BKLcusQ3gasFVtPRtCbm0bm5nf137C0sTlgIAzsw8U8Vslod2PcRvh3qGIk7TGnH5BUBA0+q/RBYIgiAIAigvsB1rQBYIrUGLpKIkmzmT/pyEFadX1NkaEvMSa3R+qFcoIDP1w2hi3UtJQBAEQRBAeb7tmJsFRFllmWB/ydElovO++u+rOluDXsy1YweD0YBKY6VgLNQzFPAwuS4KU2pzaW6HXBgEQRCEHQtEw3FhVEWpvhReCq9aX4OzAsJgNOChXQ/hfO55wXi4dzgQ0ROABMhPAkqy2EyXJgBZIAiCIAhbC0SP6YCne4P+xGIeAGBMzBibsboKsLS2KNjjcPphHM04yqebcvQI6QGoNUBAK3Ygu2YukYYECQiCIAhCKCCCY4GJy923FhP2LBBilob8ChEXTC3grAUio1S8amYH/w7shn8M+7Og8aSiVgUJCIIgCEIoIAY96b51WGAdA8HhIbdt8lVXAsLSAsE4KCVZqi8VHZdJZeyGX0v2Z9510XmNERIQBEEQzZ2yPODSNnb71ueAbvfVydsYGaNNdUZHFOuKRcfVIk2+6sqFobfInNh4ZaPdeaWVtgLi59t/Nu+EdmZ/XtzSZGpak4AgCIJo7iSsMm/HDAYkddP16ZUDr2DYr8NwJP2Iw6d5DnsCwkvhhWXDlmFAiwHo16IfAKBQW1ira+WwdGG8Hv+63XlleqG1pJWmFToHdjYPdJ0KQAJkXwBKnRdRDRkSEARBEM0dyyJMddiBc/O1zQCAD499iLG/j8Ubh95wOL9IVwQAWDZsmWA82CMYI6JH4OtRX6OVLxucWKwXFxs1xZ4bxRprF4a/yl84wcMPgEk0fTuy5gtrAJCAIAiCaO5whY4AILxHnb9dYn4iUktSseHSBnz939eic4yMkc9oiAuKQ6A6kD8W5hXGb/so2VoVJWKVNGsBZy0blgLCV+mLRf0W2U7iXC/5SbWwMvdDAoIgCKK5ozPd/HpMB2T1Wx7o85Ofi46X6ctgZNgmVD5KH0HgpKWA8FZ6A4BN+mRtUKovRZG2SDBmMIp31ORcGK8PeB0H7j2A2IBY20kTPmN/isRwNEaokBRBEERzh/Pf10EhpurCxT8opAqoZCpzNgNM1R1NeCtMAqKWLRCfnfgMK8+stBkv0BYg0MNsDckszcT1out8EKWXwgsSezEkEb3Yn7Km0ZWTBARBEERzhyvYpLBNj6yXtzfqoZAqBGNc/IOP0gcSiYS3RgDCLAxeQNSyBUJMPHDrshQQY34bg0rGnOrpq3RQvZP7/eqdi6to6JALgyAIornDuTCUdWeBsGf6B4Cfzv9kM8ZZILgbsr2sDc6FYS9jo7bhhA2HpXgAAI1KY/9kTkAY9U2isRYJCIIgiOYOJyDqMAPDUV+LG8W21RmtBUSHgA6i5/JBlHUQA8HRMaAjv20dE2GNQwuEpUBrAlYIEhAEQRDNmTMbgLMb2O06cmEU64rx8oGXbcYX9loIAMguy7Y5dqOIFRWcQFjUdxFGx4zGD2N+EMzjXBj2KkHWBne1vwv9wvoJ1mUPhxYImRKQmG67dvp8NCbqVUAsWbIEffr0gY+PD0JCQjBp0iQkJgobi1RUVGD+/PkIDAyEt7c3pkyZgszMzPpcJkEQRPPhtznm7TpyYTy37zn8k/KPYMxL4YXWmtYAgKyyLMExhmHwccLHAMyFnEK9QvHRkI/QM7SnYC4nIBy5MLLLspFS5Hwrba1BK9hXy9W4lH8JAPD+sfcdnsutRxSJxByoqqs7wVNf1KuA2LdvH+bPn4/Dhw9j165d0Ov1GDVqFEpLzb/Ip59+Gps3b8b69euxb98+pKWl4c4776zPZRIEQTRP1H51ctmDaQdtxkI8QxDiyba1thYQBdoCfvti3kWH1+ZiIPRGvc2NH2DrSQxfPxzjNo5zOk7Cep5UIkXnoM428/QicQyW2SKi8IGUZIFwie3bt2PWrFno3LkzunXrhlWrViE5ORkJCQkAgMLCQnz77bf45JNPMHz4cPTq1Qvff/894uPjcfjw4fpcKkEQRNOiLA/YMAe48rf9OVzHSCcp0hVhzo45WHdxnd05Yu4JgBUQnLnf+oadVprGb/cO7e1wDV4KL0ggEb0OAFwtuGpeS7n4WqyxDpT0kHvgiR5PAABkErNAKDcIRcDL/WzdNDY0oUwMt8ZAFBayFb4CAtie8wkJCdDr9Rg50lzmMzY2FtHR0Th06JDoNbRaLYqKigQvgiAIwooDn7CxDmum2J/j39KlS645vwZHM47i3SPvih6PT4vH8PXDRY+FeITw8Q06o05gPUgrMQuIl/q95HANUomUb+8tFgfxX/Z//LbRaLQ5LoalEOkX1g+3RtyKUC+29oSBMfAZJeUWVoSX+r2Ee2PvrfrinJuIBET1MRqNeOqppzBo0CB06dIFAJCRkQGlUgk/Pz/B3NDQUGRkiPdaX7JkCTQaDf+Kioqq66UTBEE0PkqcePp2MYgytzyX345Pjcevib8iozSDT7l05H7wU/vxN34AWHfBbMU4nnEcAHBnuzv5G7cjuOuIFZO6Xmhun60z6qq8FmDOtOgY0BHfjP4GCplCUAmzwlAh+Okp98R9sU52MCUXRs2ZP38+zp49i59//rnqyQ5YtGgRCgsL+VdKivOBMgRBEM0GryDzdpmp9XV5vnnsqbMuX9LSJfDw3w/jrcNv4bYNt+H3y78DEI8R4PCUe0IqMd+CPk74GHqjHhmlGVh7cS0AYHDEYKfWwVkyrBtq6Y16/HDenLWhMzgnIDgLBHddAFDL1LyrhEtJrag0CQhX0l+5uU0giNItlSgXLFiAv/76C/v370dkZCQ/HhYWBp1Oh4KCAoEVIjMzE2FhYSJXAlQqFVSqplEWlCAIos4wWhQ8KskEPAOAm2z8GQJaA36uW2/t1V5YfGgxugV3c/jEr5QpbcZSilPw/L7n+f0oH+fWZK+c9dkcoSiybM3tCOsaFAAgkUiglqtRXlnOui48zEJCLXOhtwUnIMgC4RoMw2DBggXYuHEj/vnnH7Rq1UpwvFevXlAoFNi9ezc/lpiYiOTkZAwYMKA+l0oQBNG0KLCwzmpNT+pcQGVU/2pdssyBH//v5L8dWiCsS1cD7BN9Yr45tT/YM9ipddirRplZKiwB4KwFolDHxudZWiAA1moCmFt8Z5ax17csbV0lTSiIsl4tEPPnz8fatWvx559/wsfHh49r0Gg08PDwgEajwZw5c7Bw4UIEBATA19cXjz/+OAYMGID+/av3B04QBNHsKc0FEreY9ytMwebX97M/Y8dV67LcjVSM9NJ0QdyANZyAiPCOQGpJKgDgnr/uEczxU/k5tQ4fBXujtw6i5G7wHFUJiDJ9Gc7mnOU7hHLChIP7PJzl4WbxTQBApE8knIaCKKvHl19+icLCQgwdOhQtWrTgX7/88gs/Z+nSpbj99tsxZcoU3HrrrQgLC8Pvv/9en8skCIJoWmRaxTdw5ZiL2BsggtpX67JiFohhUcMAAEmFSTY37PGtx/PbnAvjx7E/2r2+ZYyEI7xMN2XrGAjrFNKqgigXxy/GnJ3mwlqWJawBwMNkPeCE080Sk4DwdkFANKEgynq1QNhrhmKJWq3G8uXLsXz58npYEUEQRDOgokC4ry0CdGVABWuqh494jFlViPW3mNFpBvak7EFeRR4iDeYb6667diHMKwwHUw+iQFuAAeGsWzrYMxhxQXH4L+c/wXUe6/aY0+vgLBDWMRD52nzBflUWiG1J2wT7I1uOFOxbWyC4bJH2/i4IsCYUREm9MAiCIJo6ZbnC/YoioMSUGq/wBFQOGkCZYBgGbxx6A71+7IX41Hj2slYujKVDlyJQzcYD5Fbk8jEQz/d5HmFerEjZeudWbLtzmyBAUiW3DYQf19p5twrnalh9fjXO5Zzjx7ksCQ5ngyg5rIMjLQVEfkU+koqSAAADwwc6f1HrIEonHqwbKiQgCIIgmjrWAkJbDBSls9s+LdgeDVVwLOMYNlzaAJ1Rh9cPvQ69QY9Ko7CVdZ+wPnxAYbGuGKWV7FO2SmYWCD5KH5uYAcvjlvOcxbL/xLyd8/htrk4Dx5env3RaRKhlakisfi+WAoKrgaFRaWxiJRxiGUT56wPAF/2BStsS3I0Bt6RxEgRBEPVEdiLwz9vCsfwkoNhCQDiBZXZEeWW5wPrwav9XoZAqoFFpYGSMkEvkqGQq+SwIsYwLSyzLQ3NolA66WlphKTYs4yC0VjfmjNIM/JP8D0bHjBa9ToA6AHkVbI0Ma/EBWAgIfTk/j7O4OA1XjyP/BpDMWnKQmgC0dMGK0UAgAUEQBNFUKc0RdtvsMw84thI486s5kNKJ+AeGYXCt8Bq/7yX3wurzqwEASqkSd3e4mz8mlUihUWmQW5HLN8kSq/lgiXUsxYYJG6puSmWBZUVLwXUNtjEaueW5MBgNOJ55HJ0DOwusB9YWFWu4NM5jGcfwwzm2QFWAOsDpdQIAIvuyPznxADRaNwa5MAiCIJoq6+4DMs6w212mACMXm49d2s7+rEJA/HDuB8StjsOGSxv4sbTSNHz939cAxDMbuBs611WzKgFhWRL7z0l/okNAB4fzrbF2d3C9KqwtEABb/GrT1U2Yu3Mu5u+ez48zDGO3MBYHZ4HYe3MvsspZceRSDQgACGpnbunNsWoccPO4a9dpAJCAIAiCaKrcPGrejh4AqER89b7hoqcWVehRaTDio+Mfufy2XGlnBuyTtVLqWEDkVOTw2601rV1+P8sYCAB8cKOYG6JEV4K/rv0FADiRdYIfL6ssg5Fx3GxLrK6FZbVKp5BIxLuebn/Rtes0AEhAEARBNAc4S8PUH4TjGtty0dnFWsQt3ompX4l3Qa4KztTPoZA5joEI9nCu4qQ9rIMYL+dfBmDOwnihzwv8sRJ9iSADhEv9FGsFbo2YgBALAK0SMQGhtxU7DR0SEARBEA2EgjIdbuQK6wOcTyvCnFXHcD6tqGYXD2W7HqPTRMDypifSA2P3hUwADE5nJNocsywGZQ/r5lLWgsKa9299H7dE3IJ149c5nGcPawsEV/+Bs0AMCB+A5/uwPTb239wvCNrMrWDdJ5fyL1X5PrUmIILa2Y41wsqUJCAIgiAaCH3f3Y0hH+5FaoE5+G/Gt0ew+2IWHv0pwfUL+pjcE8NfBQJMvYckEraRFodfS5vTPJQyyH3OwqvNUptj0ztOF+zr8gaiTCcMPrQWDF2DujpcZnv/9vhy5JfoEtTF4Tx7WFsguNgLLgZCLVfzN/rMskxcKbjCz82vyMe5nHOCeAgAmNx2ss371JqAiOhpO1ZJFgiCIAiimugqWR/88aQ8fiy3lA1SvJFbjSdUrtphp0nCcVOQIQDAw9/mNLlUCmXwLpvxuOA4QYyCoTwC2sxxuJYttJpYWiDu6XCPSxkV1UEhVWBo1FB+v6CiAAzD8BYIlUwFX4tiWZcLLvPb+RX52HhlI7/fJbALdt21C4sHLrZ5HzEBUVWAqCghnW3HSEAQBEEQ1cFgNKfyMQxwIjkf0785wo/JpFUXexLAMOZUTZVVUSbLYkoiRaSKteWQqbLMUyDFyRkn8ePYHwXiwFARCUCO8+lC94plOqRLfSJqwOfDP8ezvZ8FwFogssvNfTA85Z4YGW0uS20Z75CvzUdSYRK/7630RphXmGgfjlqzQIgFrjbCVE4SEARBEA2AUgs3gJFh8NnuyzhwxZydoJBJsOW/dGw9k+7cBbcsBExZEDYCwuC43kFSsTD2wUceALlUbnNTlUjY6/xxMlUwblntcVTMKOfWWwtw3TsLtAWIT2PrLHQJ7AJPhSfkUjna+rW1OSchMwFHMsxCzVEFTK6ZliViZbirRCkWE8KwLdcv21p+GiokIAiCIBoApVrzTb1Cb0RyntBlUaE3Yv7aE3jspxM4aCEsRGEY4Ph35n3rG98dn7I/h4inDhZohRaFYEVH0XlGbQgACGI2ALYRVr+wflg9djXCvcXTROuCYE82m+NKwRW8evBVAEDXYHP8hZi7YdPVTYJ9R2mZtWaBEKM8H1jWBfjpLrYyZSOABARBEEQDwFJA5JfpkJJnP+bhTGqh44tZ9r4I6mDrpugyBVh4ERgqLiCKdUIB4QlzoGVhmR4vdv8c2pwh0OWz5ZfzSoXFpFr7tcY3o79Bj5AejtdZy3DuEq4CJgCEe5kFTFUltQH7VS2BWoyBAIBJX9o/lnu1etesZ6iUNUEQRAOgRGsObDydUgC9wb5PvExnsHsMAFCQbN6e/pv4HF/7PTBK9EIB4WWMBcBWa+z25k7T6Fgo5VLoKo0orqiE3mCEQubeZ9IWXi0ggYQvYAUIXQzO3OwdWRRELRDSalogut8P/P2GuSuqJSVZtmMNELJAEARBNAAsLRCHr+U6mAmUaR3HMODGQfZnZB/ROg9VrqWSLa5k1GtQnnoPJLoIpOSV4fN/rgjmhfmqeeNGfpltSevqwDAMvtx7FfsvZVc92QqFTIFWmlaCsS6B5tTQqipiAmzKpz3E6llUKwaCQ2untkdJZvWvWY+QBYIgCKIBUGIhCooq2O37+kZjYvdwfLgjEQk38vnjZXoHFoiiNGDnK+x2zC3VWguXpaAv7I7Koh7YeT4Tx2/k27gqvFRy+HkokF+mR36pHiE+9m++zrLlTDre334RAJD0XtVFq6wZHDGYb/wVqA4UxEBYVsSUSqSipatdtkDUJAbCXvEoskAQBEEQVZGcW4biCr3AAsHRsYUP+rcORLdIP8F4uSMXxhlz0yv0nFmtNRVxMRBG8w3TWjwAgIdCCl8P9qZcXKG3OV4dDl11bH2pihbeZtfM2FZjBccsLRAB6gDRm/+A8AF2ry1mnaiRgAi3EyNSkAwYHfflaAiQgCAIgnATSTmluPXDPRjx8T6bYkwAEBvGZgS0DxVWWhQTGzw5phTMYS+bq086ycbLG7H7xm6UmbpS3t7FcWMrT6UcHgq2SFSVcRlOkpJv24LbFbhUTsC2QqVlDISP0gdag7lb5667duHX239Fe//2dq8tlUht2ndbl9F2iak/AH0ftg2oTI4HPmgF7FlS/WvXAyQgCIIg3MTfF1hfd1axFv/bc8XmeIdQtibBgDbCltHljlwYpaYneC/XGlSllqTitfjX8NTep2D0ZFuAR/gGODwn2EcFT6Ws6jU5wfaz6Yi/koOKGl5HICAU9gWEdYGrMK8wdAwUT1e15JOhn6BPWB9+P8gjqJorBeDfEhj3ARAmUuq7ogDY916DLjBFAoIgCMJNVHVv0Hiy7oGWgV5YNDYWagX7le3QAlFqCj50UkD8EJ+Ej3YkIrvUNmixd7TjGg4hvip4cAKiBhaI5NwyPLLmBO7/5ggsC27qDa6b8S0FhHWXT8s0zpa+5tTUfi36OX39XqG9MKCF2c3hKO3TaQSWEquU2xOra379OoIEBEEQhJtIK7RvrrcuXf3wkDZY+UBvAFW4C8pMRaa8qn4yNhgZvL7pHP635wq2nku2OR7o4Y9/nhli9/xQHzU8FPKq11QFlzLNpaUtKnpXyxph2fOid1hvwTFLAdHarzWWDV2GMTFj8NGtH7n0HoVacx0OiUgpcJex7F5q3Whr8xM1v34dQVkYBEEQbiKrWGv3WKiPbXCej5oLWBSxQFTqgDV3AvlJ7L5n1QKiqNwc+PjnmcuAn9X7KX3Q0tcb0QGeNpUxASDUV827MF7aeAbDYoPRQmObqVAVeRYpoFxDMYCtvulqYkekdyQmtJ6AQI9AhHiGCI6VV5oF2y3ht6CFdwuMaDnC5fV2C+kGnHf5NPv4hAK95wAyJWCstK1EaTQC0ob3vE8CgiAIwk0UltlmLoT6qpBbosOye20j9AO9WB9+dokWDMOYn35vxAN73wOS/mX3NVFOBVAWWAiIvIp8WN+rubLO/l5KUQER4qsS3PD/PJWGR4a0qfJ9rbHM8LDM5qiOBUIikeDdwe+KHksqSuK3LbM1XGVk9EgsHboUnQI7VfsaNtz+CfvzyNe2x/54BLhTZNzNkIAgCIJwEwXltqmRz4+Oxfi4FlArbFtgB5usErpKI4q1lfA1WSTwvTBdET1mAE600C60EBASmVAgGEu68fEEGg/xEtChPmpBAan3tl2EWi7FrEGuZX8UWAipIgvriraydjI7OOZ1nYcz2WfwQt8XanQdiUSCkS1HVj2xOgS1sx3775cGKSAank2EIAiimZBfyt4424aYg+jahXqLigcAUCtk8Faxz33ZDtwfGLjAqfcvsLj5S2TmNFKJPhQrRi3lLRweCvFbRYivSnDzB4DFm89DV2nEoau5TlsQLIWCpVulpjUhrBkUMQhHph3B3R3urtXr1iotuomPV9FB1R3Uu4DYv38/JkyYgPDwcEgkEvzxxx+C4wzD4LXXXkOLFi3g4eGBkSNH4vLly/W9TIIgiDqHswC8P6Ur1AopgrxV6BKucXhOkLfJjcEJCOuCQ1H9AKVzmQH2LBBv3focBrczZzBYipXHhppdFGqFDHqRgkcTPj+A+1Yexqt/nHUqO0Nr4Qax3H71z3NOfArXkEsbuOHdMwC4ezXQ8Q624RlHUar9c9xEvQuI0tJSdOvWDcuXLxc9/sEHH+Czzz7DihUrcOTIEXh5eWH06NGoqKio55USBEHUHXqDkS9f3TrIG38vHIIdTw2GVOo4qj8qgI3YT8oxWQys+ykMf9XpNfACQlIJheYUe7nskRjf5jbBvJ7R/gCAEB8Vnh3VAff2icLbk9geE+/dGWdz3URTVsX6hJu488t4GI2O81W1+oZfdbFe6TQRuOdHtuFZUAd27Oo/7l2TCPUuIMaOHYu3334bkydPtjnGMAyWLVuGV155BRMnTkRcXBxWr16NtLQ0G0sFQRBEY8by6d/XQ4FIf08EelddFrlNMOvuuJrNVotEhSmlUKYCXs0FWg12fg0m94PC9yQ/1i8iDnKrrppPjmyHF8bEYuuTrMB5b0ocpvdn6yj0bRWAr2f0svseF9KL8Pomx5YER7EOJ5Pz7R5rFnQ23Sv/egr4tDtQ4nqTsbqiQcVAXL9+HRkZGRg50hycotFo0K9fPxw6dEj0HK1Wi6KiIsGLIAiiocPFDviq5TY1HxwRG+aDYdKT0J/ewMYYcALCww+QuWae57IwJIoCfuzhwbE283zUCjw6tA2C7Agce0GWHD8evuHwuKXbwprJX8Q7PLfJM+hJINRUqTL/OvD3Yrcux5IGJSAyMti+6KGhoYLx0NBQ/pg1S5YsgUaj4V9RUa63riUIgqhvuABGf6+qW0xbMrZTML5XfojFuo/x1pc/oKTQFGio9nN5DRmFrGuYMZiDODsG2u8FYY9Ab9c+gzWOBATAWqebLUpP1p0R2Jbdzzzj3vVY0KAERHVYtGgRCgsL+VdKSoq7l0QQBFElnAXCr4qnd2s0xgJ+OyTrX3y34wi74xkofoIdsooqsOVMOrsjYddiKGuJQA/XrgPArmXCWbRVZGvUVqOuRktAK2DqKnY7/XSD6dTZoAREWFgYACAzM1MwnpmZyR+zRqVSwdfXV/AiCIJo6HDuA42ni0/vpVn85pPy3xGYfRQAYPSNcOkyey+ZfekSqUlAaEPtTXdIVS4MQJgyak1VFogSR70/mgsai+Zfu99w3zosaFAColWrVggLC8Pu3bv5saKiIhw5cgQDBtjv0U4QBNHY4G6oztx8BZzZINidJme/L88WeYrNtkuZ5U3ZZIEA4+JauNMt+kHEhvmgd0t/mzmr4pPsnl+VgBAt3d3csHRRHVzG/vxxMhtYeUM8RrCuqXcBUVJSglOnTuHUqVMA2MDJU6dOITk5GRKJBE899RTefvttbNq0CWfOnMEDDzyA8PBwTJo0qb6XShAEUWfkmso3B7oYA4H4z0SHTxZU3YPCYGT4ugxppviH1sFe8DOVjWCM1RMQgNmN8dakLvhwqm0xpGV/X8Zli6ZZllhnYajkwlsTWSAASCTAAoseGRWFQP4NNrDSTdR7RY3jx49j2LBh/P7ChQsBADNnzsSqVavw/PPPo7S0FA899BAKCgpwyy23YPv27VCrXeyoQhAE0YDJMRVnChZpmmWXcvspjVqln8NTP999GR/vugQA+OeZIXwA5f19o3GsxAvxWUDH0ADn12LFrqdvRXJeGbpF+YFhGMzo39Im+2JVfBLemdxVMKarNOJadqlgzNoiUUIWCJagtoBHAFCeBxSkADpTKq/K2/F5dUS9C4ihQ4c6jKiVSCR488038eabb9bjqgiCIOqXnBJWQAS5ksGw8VG7h1LLhF/nDMPglT/OwketwAtjOvDiAQBW/nud72Hh76lEACRAFnBHtxjn12KFv5eSzyiRSCR4a1IX6CqN+OW4ObBdJbct0f3Xf2mi1+saocGZVDZFdfq3R3D5nbFQyBqU1909aCJZAVF4E9CaLDoqH7cshf41CIIg6oCsYsfVc3NK2Bu4SxkMDszVScUSQcXHxMxi/HQkGSv2XcW5tCIABsh9T0IiL8C6o8m4kM7efDQeclwtuAoA8FK43orbEe/fFYcP7jJXqlTKbW85lgW1LPl+dh/B/vGkZl5QisMvmv1ZcAPQm8qPK0lAEARBNAlWHbyOvu/sxho7BZQYhkFqQTkAIMTHSfesQQ/ksjd6jHrH5nBepQqZFqLlcmYJv33nl/FQ+B2DR8Qv8Gq9DIDZAvJv9s+4kHcBAKCS1SwdUxQLg7OY9dlelesgbxXGdjFn36nsNPRqdnDZGFkXzGNucmHQvwhBEEQts3jzeQDAK3+cFT1+M78ceaU6KGQStAt18su/NBsw6gGJDPANtzlcAg/czC/n969kmQWErtIImRcrPiQyoWXk96Rv+W2PWrZAAEBclLk5WJFVLAPDMPj8H/vNEsfHteC3DVX002gO6A1GZEiC2J2E7wEABokcGaXu+d2QgCAIgqhDxJ6695lqMHRq4Wu3dTfPrteBX2YAZaaKkx7+EDzWmyhh1MgrNddasNwGAMZgmeYpnjYZ7mUrTGpKbJgvxnVlLQlFVu6KQ9dybdqBA2xmCADcHmdej66KVM/mwCM/JuD1/SWCMRlTae6LUs+QgCAIgqhFKqyqKlpnImw9k85bJrpF+VV9wYPLgAubgBW3sPse/ubeCBaUwkMgGvLLrAWE2VUikbPxD94qYeBlS9+WVa+nGoyIZQtUFVUIxcLFdPG0zp/m9uO3O4ezxQF1BtcExPGkPEEb8saOwchg98Us7DN2Q65cWPDLR+2eFuUkIAiCIGoR67iHTyyyHwBgX6K5AuSgtkGOLyaWsebhDwS3B2ZvAyZ/zQ+XQcULCIORwV//sWWqQ0xpooEWcXYSGevqGNVJeCPSqDSoC7xMQuXfyzk4lVLAj1sGULYMNFtIWmjMrhQu8FLvggXiwOUc3LXiEG7//N/qLrnBUWwSXxVQ4b3wz4FBT0LHyPFj5Uj4qqtfv6MmuEe2EARB2CP1BHuTDGjl2nkVRYBE6raAMo6TFjdIwLbSZEo+Gzl/X99omxu4DQaRDAUuZa/lQCCaAW4cxIEMGXBNgl+OpWD+sLZYb5E6+cFdcRjaIQQvH4jHJlMMpkTKxkEEaEoAU0jEd6O/c+4DVgMvldlNM2n5QSS9Nx4AkFZgjtkY0j4YA1oHolO4sB0Bl7rpigVi48lUAEBmUdOxQFi6etKNfqgcvhjtd/cFIMEJV6uZ1hJkgSAIouFw4S9g5TDgy0FA1kXzuK4USP9PvInQhc3ADxOA96KA5X2B4kzbOfXIzTxWINzZk+1NYR2LkGxx3LIEtCgGkRugxOJrWyIB7vgMlzs/wV87u1iLSxYZGJ5K9jlx09VN5vOk7HWLpKcBAH3C+qBPmDBtsjbh1mDNjdwyfltvMGJs1xZoGeglmKM0CQi9CwIit7TpCAeOAgtrTX6ZzlSdk/37IRcGQRDEuY3sT30psP1FIPMcu//LDOCrwcCmBWzd/3X3Awk/AL/OBH6ZDlzfz84rSgVOr6vzZeaX6vDkzydx9HoeACAxoxhJOWw1RS49c0pPNt2uuKISeoMRFXoD8kt1/FN3dIATvSsqRRpQDXrSZmhKL3OjpfirOcizuIH2aumPK/lXBPPv7hOM2+NaIB+sgBgSOaTqtdQASwsEYA4sTco1V6DUVYpnEphdGM5nGuSW2G/c1VixdPcUlOlRVM5mtHgqZW4rsEUuDIIgGg6ZFmmP1/YAXw4EZCrzk/ipn9gXACRuEb9G4lb2JlvV030NePH3/7DjXCb+PJWG/c8Nw4T/HYDeYMTk7hF8gajYMB9IJWydg7xSHaZ8Gc+nWXopZXxsgkOsLRBjPwBaDbaZ5qtWYP6wNli+5yp+OpLMl4JeMb0nZFIJcityBfN7t/bElDE9Mea3ZABAt2Db3hW1iZeVBUJnMKJCb0SWRZDj/GFtRM9VyCT8Oc6SauEaaSpYdjPNLtEiz7TvrvgHgCwQBEE0FPKuA9mJ7LZlZT0xM741t70JdJzAbqccAd6LBpIO1v4aTRy4nMNv/3g4CbpKIxgG+N3keweAAC8l/Eytuo8l5QlqNLQO9q7afQEAlVbVLL1D7E6d0T8GUglw9HoeTpviMFoFsfEgJTphml+pvhQMwyC3nBUWQR5VBHPWEE+l0AIxeul+fPvvNQBAhJ8HLr8zFq2DxWNX+BgIJ4MojUZG4DZasvUCtp5Jr86yGxSWDcV0lUZsO8t+pgBXm7HVImSBIJocxbpivHbwNQyPHo4JbSa4ezmEs/z9OgAGiOrPFko697vweL9HgfDuQFke4B/DFtLpMA7ociegNmUPxH8O7HwF0BYBfzwCPHEKkDqus8AwjHM3cxOVBiNKdeZUzZX/2paXHtslDBKJBCE+bGbEjnPCuIxWQV4254i/mZUpnitjLEKYRo0e0f5IuMGWfJZIzJkNxXphumSpvhSl+lJUGFiBEugR6Nx6qomXVbpoUm4ZPvuHdat0CPNxaILnXRhOWiC4GBOOr/azQoUL3GyslGmF6cHfH0gCAPRvXbf/do4gAUE0Gop1xTiecRy3RN4ChdS+2W7lfyvxd/Lf+Dv5b9ze+naXbg6EmyhKB87/yW4PehII6cg+fbccBOx5B4gZDAx6QliBMXac7XUGPs7O/WYkUJAM/L2YtU6I/A2Uaivx5ubz+PtCJhbf0Rnju7aAVFr138rFDPHaBfzxt8bw7ajD/TxwMaMYh68JXQhcoaQqsba+aOwLCABoH+rNC4hukX58kSprC0SBtoB3a3gpvOAhr/0KlJZYt+e2xLeKAEBXgiiziisw9KO9oscKy/U2GTGu4qrYrE3KdEIBwbl02oa4L+uIXBhEo8DIGDF7+2w8secJ/Jr4q8O553PP89txq+Nwo0i8HwHRgNi0wLzdegibwnnfOmDgAmDRTWDar6Llm0UJ7w6MXMxux38GvBkInPxJMCUxoxgD3/sHvxxPQW6pDo+vO4kv9l4RNKMSo0JvwO2fHwDApmd6WZnmA7yUUCtk/E0m3I8t3mRd0Mieud4GawuEl2NXQxuL63JZIAArvgFALmFv1tcKryGnnHXDBKrr/gnW0U1XrEOnJeY0zqqDKN/bdtHuses5pXaPOYJhGPx05AYWbzqHrot34lhSXrWuU1PKdOItza0DVOsTEhBEjckqy3LYor022JOyB4n5rH98T/IewbF9KfuQUmzOe08qShIc//DYh3W6NqKG/PM2cOVv877S6um8CheEKAPmAx3vYLcZA/DnY8B/6wGwImD0sv02XSA/2nkJfd/djYNXcqyvxrP7Qha/HRXggX+eHYozi0ch6b3xWDOnH7Y8cYtgfudw8cJM/VoFOPc5LC0Qs7dVGRhq+X6cmySjNANfnP4CANAjtAcA4HD6YcSnxQOoe/dFVVTVJItzYTgTA5HowDqUWeS4O6o9Np5Mxcsbz2JVfBJKtJX4at/Val2npnAWCGvRah2gWp+QgCBqxOarmzFi/Qh8fvLzOnsPvVGPp/Y8xe8fyzyGpQlLcc9f9+CPK39gwT8LMO73cVgcvxjfnf0OmWWsv3l+9/kAgH039/HtiokGhrYE2G8WeBXjP6ud60okbHyEJXveBhgGvxxLET8HbIfKad8cQcIN26fMSoMRy/42V5Us1RoQ6quGjykK/pZ2QYIKigAwtVck/D1tzeahvlV04Pz3Y+Db0eb+F6Fd2cJRVdAj2o/f5gTE9uvb+bGhkUP57a//Y6tY1nUAJUekv7ibxJF7AzBbIKpyYRSW601ty8Upt3IBOMs/F7ME+yeTC9xSIpsTEFFW6b+eZIEgGiNFuiK8dOAlAMDKMyux+ermWn+PisoKvHXoLcGYkTHiu7Pf4Xzuebx68FV+/LfLv2FpwlIAQCtNKzzS7RH0Cu0FAHy74qZAbnmuwOLCMAxOZJ5Amb7MwVkNlOv7BLtf36jFZk5B7YT7+UlASSbO27nJtLPwJR9PysfexCx8sjMRp1IKsGLfVaxPuInLFh0uh7QPrnIJcpkUuxYOwaKxsZjePxptQ7yx4ZEBVa9995tAymHg0BemCzkXaa9WyLB2bj98Oa0nIv09sSd5Dz5O+BgA+3/i/o7325wToHbSGlJD1s7tLzpelQtDaUrjrEpAVNVQyjqGwFmsU0JzS3WYuiK+WteqCZwLwzrmwV6RrvqAgigJpziZdRJP73kafcL64OG4h9HWvy2e2/ecYM5LB15CgbYAMzrNqNZ7xKfG49n9z+LJHk8irTQNiXmJOJF1AuWV7H/gMTFjcGvkrbxoccSAFuyXdKR3JBIyE3Ct4Fq11tTQSClOwZRNU8AwDNZPWI8YTQw+O/kZvjnzDToFdsLP439uPEGjhkoYzm0Cd/u4yQThuzM6zLpdXzu57aFd2JfaD8i9DJRkQnvyF+y60AkA8NbEzlDIpHjx9zMYHhuC72b1wZM/n8SuU9cQlfgdFm7rgjz48tkCAOCNMiyLOYKMVpMwcUh7p5YR5K3Cw0PEaxxUSZ7p71bmRM0IEwMt+ms8secJfvuONndALpWjlaYVrheaM0fqywIRHSheOMtZC0RVLoySCvEYAQ57MQRVIXbdpNz6F+ucAOoSoeH7nAC2Lo36hCwQRJXsv7kfD2x7ALkVudietB13broTp7JO8T5USz449gHu3nw3MkozqrzuD+d+wP9O/g9Ghv1iePjvh1GsK8bbR97Gd2e/w8G0g7x4AIDuId1xe+vb0TmwMwDYjRxv798ec7rOAQBEeLOBZCvPrITWmXoCDZz9N/ejvLIcFYYKbLq6CT+c+wHfnPkGABs8ej7vfBVXcD8nk/Nx79eHUPztBMjO/AwAeF72HOZ4f4ECLbDjbNV/O9YwDIPF8YuxLGGZeVChBh45AMzeAoSx3StV/7yG/NIKhGvUmNo7Cnf3jsJvjw7AVzNYS9WgtkF4Rf4jxqX9D18ql2Go9CR8wQbf3SI9g7PquRiZsRLTsz6Gj1LGVso0Vu/J1ilKTeZzJy0QjvBT+QEA/jf8f4LxRhMDUYUF4mxaocPj1XVhNBS49Uf5ewpKV3uqyAJBNGBWn18t2GfAYMY2s5Xhm1Hf4GjGUd6neiHvAp745wn8OsGcLbHu4jpkl2VjQY8FkEqkyCjNwEfHPwIAGBgD7u1wr8M1PNDpAUxuOxkSiQSrx67GobRDGBgxEEmFSfgn+R9M7zQdxbpihHqGQlKaDfz+KBB7O3pE9OCvkVKUgrb+bWv8+3Anlk+OK8+stDl+MPUgL7AaKpO/iIc/iuCjNgtQr85jMdrHB4m7L+Pn04fRrz2DaN9oMAwDbaUR13NKYTAy6BJhG5RoZIy4nH8Zv13+DQAwvdN0/qk6q1iLNzafh3f6ULwPNlBzpPQEhoX5QZ1wGfAMRK+ud/HBiZ1beKOznA3S7Se9iH7Ki9ht6IHH9E9ijXKJ+U2v/gO86c9u3/E/oGf1rG6iiAUke1XtLqkKfzW73nBvoZuoPrIwHOFsFoa+iiyMD7Ynio53i/LD6ZQClOnrT0CcTS1EVIBnjdNGLSk1WVA8VTIEe6tQbLKMkAWCaLDklufiWMYxAMDCXgvRPbi74Pjd7e9Gvxb98Ei3RwTjF/Iu4K1Db2FP8h6UV5bj3SPvYuWZlfjzCpvrfyjtED/3mzPfYOSGkfy+TCLD832e5+MXAOC5Ps/BU8GaQJUyJYZEDYFCqkA7/3Z4uNvD8FJ4IcyLLd6DHS+xUf1/PYX+O99BoJK96XDBlY2ZpMIkh8cPpB6on4UA7I3uwma2L0XaKadO4Z6ipsjMbZZ/MwzGoNhITOweDokiD5fkb2LiH5NhZIx4+pdTiH11O8Z++i9u//wA4q/aZkh8cOwD3LX5Ln5/6b87UKKtxLGkPAz5cC+2nEnHL7mtscvA/j2tVH6C+2+8Bmx/Afh9LpC4jT3RUIlO38faXH+E7CS+Vnxi/0MdsxVyNcIg0sch2HZdrhLuxQoHuVQOH4W50md9uTDsUaULw4l23pUi1olXxnfE2nn9MNjk0qltC4S9lN8j13Jx++cHMOXL2o2TKOezMOR8fQ8A8CABQTQ0SnQlmLdzHob+OhRGxojOgZ0xu8ts/DjuRzzW/TEAgEqmwjO9nwEAKKQKPN/necE1fr30K57Y8wQ2Xt7Ij/2byt44/rjyh+j7vnvLu/j33n8xo9MMPN/necilckzvOB3QVwCbnwS+GgJc3QNUatnSxwDbxtlQCTAMKsqKUXhup/mC1/cjpDgbALBg9wKRd2xYMAzDp8QaRL6g0nLOiZ4XVsk+jVzMPlN3i7OkUgu84cc2skrcAnx7m+P5ps816/ujkMKIufKtAIDTxtZ4A/MwoE0g2gR7o2cHVuRVMjpczL2KP06lCS6zIeGm6XLm381PF4Q1Hn49dQZdXt+BqSsOoVxvQLCp58QRo52b8DbT321qAiSm0tEnjUJL1RDZfwAAfZvbwHVA5Ek/bS7BXRuIBcMGd6jWpTzlrOiO8I5AbID58/uqzC2zG7oLQ+VEO2+tiLi4s2ckBrYJ4m+wpdrqxUBU2hEK2+y42n4/wZYzv5LlOKjTVXgLhFImiOeoyoJTl5ALowlwKusUrhVew4TWE1CiL4FGpYFUUjNtuOHSBhxOP8zvj4IX8HEsENwBjw58HOMn/YUw7xZQypTszV2uwoxOM+Cv9kdKUQqfdw4Iv+B33diF2dtn40TWCdH3nRAQByi8AaMRndSh+HfyNnik/Qd82Abgqun9OMl8QvQAINlszVCbXgCwSD8HL3r8icjKSlxQKVHJVEJv1DusYlkVDMPgRHIBALbLoasYjQyKtZV4f+s5pBfpMLF7BCZ2D0dReSXmrT6OE8n5gi+sIe2D0TVYikcy3wAT0hkZ+mJAIkGEvhKt9Xq8lJuHTd7emFhSgjFREShnKlF57Bv2P7bCC0g9DvSZy1Z2dH2xZoEw8AlhRceLfwnnGnTA6Z/ZMtOZ59iCR5pIVmhc3wdc3gkJgEDdE9ig3IowST7yGW8sDvoEX42P40sdR4UV45KpncQdX/8CoIfgbX4/kYp20Tn4+cZbeLrnMzh0OsZm2RKpMN//u5l9oFZI8fwPOqDMQmzc+hybQlqYwrokfryTP9R29lf47+JJxB151jw/LA6KGRvYolRHv2KFA8fyvsBixz54p9GLNILyqV52iqfCE2WVZVg6dKkguFYmMd903OHCkEsl/N95lS4MedVZGJYBlo8MaYNHh7SBxpQ+y/XhqK4Lw57lYv7aExjbZZxN9dKiCr3o/JrCBVF6KmWC1t7uhARENSjXGZBVXIFIf0/InCh9W1cYGSM+P/k5H0T3evzr/LEPh3yIMTFjqnXdzNJMPvULAMaqI3DXid/YtoLF6cC1vYgObAtoooDUBLbvQEgnIKgdbu8/H5BVoKzTTKw6/wMAILk4WXD945nHAQBBUm8MN0bjV7CBfyu94oBP4wAABqkSMqMOVdbrsxAPlqysHId1hhG4WBqNZXgLu7zYJ7G0ohS09Gvt6q+EZ82RZLz6B9sxcu4trbBwVHvIpBIwDLD5dBr2JmbjmVHtkVZQgZX/XsObEzsjOsATx5Ly0SHUBy//cQZ+51bjNfkaPKFfgKcS+2Dx5nMoKBP/Qth3KRsRV3fDW7Ef2ekHoI+OhJQBvvQcheT0bPzKhKF/xRX4TZwDnHgFAFC67TlojBZftse/Y7s4thkO+EawwYXOkHvF3PEy/jP21aIbMOBx1vRvzcaHq7zkF0pznYdTUQ9g3cxbBeZYtdJsvpcq2RoId/aMwEd3dcPQj/YiOa8MX5x/BZCV4tX4l6HL6w+lVRaiRGa+AT86tA26RPhCIpFg4/OTgYtK4GdTKuPgZ801KH6cbL5An3nwadUbca16A75lwK7X2PFbTVlHPaaxr6v/CM/TFgMqiyZg1UUnYoHwrl4MBBeEzLn/OCqN5idYtdzJv4da4OVxHfHlvqtYMb0X7v6K/b9bVXaFM1kYlhaIZ0a1F/TW4AREdV0Y5Q6Ex4r9V/HYULO1qkJvEFgmaqv0NcMwFgJCjgcGxOCz3ZcxqlNoja9dE0hAOAPDwKgrh0RbgE+PluCrfddQrjegdbAXHujfEt19ihCZ8CqYqKHY7TMFYX4eyC3RoV2oNzqE+dSqianSWInD6Yfx6N+POpz33L7ncCz9GGZ3ehQhXgF8JDPDMNiTvAfZuSG4ninHyeR8DG4XhNvjwnE9txhnS37H12e/5K8zO7U1Fur22r5B7hX2xZF1nn2Z+hlMlvphVUuzmdRf5om3/XvjpaJTKNQVQcJI8UJGEkaXnscYtQrtdXpojGahITMK/cAVjAIzdS/iKNMB82V/4i7ZfiQyUTBAipHSBNzwH4DLuToUM55oGxuHeyc9j3fe3o+TTDusKb8bHbT/IFGlRNLRL9By1EfO/roFMAyD7w6YgxjXHLgI6aHP0FKShVWG0bjMRAIALqQX4ZqpbO6QD/diWr9o/HQkGWG+amQUVeCyajUUEgO+Vi5FTMVaeJWl4Un5VhQw3ogavQCXSjygNxjhoZAhv6QC95/ZhZeDAqAyme2VlWoMuziKX8dnGIE7zoZDJVVAa9SjxMMXmtICi4Ubga3skzTjEYC88SsR2MUcc2KXjP9sx9JPC8XD/b8CkX2A78cC2ReBoPZswaPCVKAsB/AMRJFve+w/loDbZUf4084Fjsaw2W8CMuH/jVK92ewrVeThg7viMLlHBKRSCTY/fgsmfH4A+TJzSWJlwGFYE+oHbHtstE0DJwBA7Hhgzt+AZwArpEa8xtZcsMTSXdDtPiD+f0B4D6DTHcJ50VaFnVKOAG2d+L1WhZgLw8t+F057MAxjFhBy+wKiPpl3a2vMHdxKcFN11EgLMPfCcOzCMFdptL4eVyehummc1sIjxEfFtyHfkHBTICBWH0qyWpdRIJCri85g5F2anioZ5g9rg26RGvRzYyMtoAELiOXLl+PDDz9ERkYGunXrhs8//xx9+/Z1y1re2/I0/s1gq7kZIUVItApyGCCHAT+el+NniQ6pCjk8Uq4iumI1EouHosKoASQGeHtnIcDnHDwUHgjyCEWAJhyDQsajZ1g3qFTlaOHVAoXaQhgYg40v0mhkePPYP8n/4NdLv+JI2hFUMsL/COPLdHghOxtvBfnzT9oAG4Pw66VfUZ56LyqLukMll6JdTApuKD4HU+mJshuPwKgLwbH00/js5HkoAw5AIjU/Cd9TVMyLh12GnnhYvxASMHi0bQEejS2Hl1IKeAay1oeEVcARs/BobSzAj2llmBEeBgD4X8o1xF25iF88OuDv4g64V7ELKpOlvk+FOb0y2RiMZCYEPxuG4xTTFo/INuGC31D8lGPOo/+fYTJSus7H6ZQCcz62SfQ/PbI97hnJFhD6akYvPPxjAlYaxiNOdw5Q5SPp7M9QnU9HebsJSJB1xy8JN9GvVQDmDW6N6zmlGNExhK8saM1Tv5zC9ZxSeChkmDGgJXoeehxjZGyA6TT5bqQxAZihW4SrORGC8346wgojQ1EGXpf/CYXE/IWUpLYq7JNwFHhoL+DhD/zxGHBlG7Z667HJxxzoVqY330w6h/viXFoRNp1OQ1hXL2iNBTg8ZgNifNsgOkCNHG0SOpzfBmnCKqDoJiTlefBefw9Sj45DRFRrwKBHcZcHMGdTFobEhmH+cFNtA10pcLSK4MAudwFtRgAyOfDIQaCiQLRXw72f/ovz+luQw6zCLPlOPKN7BPdOeAGQ2f6ei3TmIk+xUXrc3TuK39d4KPDgoBh8fMXmNAHdW6rExQNHVB/ztlhwomXHS+8QYOEFQMwlqFADD+4AvhvN7m95Brh/PRDsXH0Iu4i5MJy1GlmgNWj5FGnrlGfr75D6hBMP38/ug0NXczGqs+OnaC6I8mRyAdILy22qfQJAhZ79nCqRm3VNLBCVBqONcPn03h64byUrXC3bhgPAuqPCKqdFFfpaERCWnTg9FTLIZVKM6Ohe6wPQQAXEL7/8goULF2LFihXo168fli1bhtGjRyMxMREhIa4r8ZqSeyMeyd6WX3bcfz4pACO4X2O5VIpETwPguZv3w1cCyGIA6IpwQ5cJFP6HXcnbYY2EkeDe8OX4+4wWShnbkOdsaiFaBnqiBNdQFrjUJnbr1rJyPFBYhF4VWsgBfJJljlDf7O2Jl4LZL3OPiJ+h974IfVEcLpdcgNIfkMjL4NVGPLL8jnwJJlVkoJfpxn7WGIMv1HNg1LP/kf93JRCJylCsfKA3AGDPxSwsPjsat/edjvSz++BdmIg3FT+gjc4sRrjtiPJEzJQnAgxQrg7B+ZbT0CuRrR75Y8t3sIvpi+dHd8CRVceQXazFK5VzAKvAe0+lDC+P64gQXzV+PZaC539jn5T7tQrA48PNTwOjO4fhrUld8OofZ3FR2x0qnz1IUsgxM/cv4NhfuFI5DnmV9+Gv/7R8YZZ+rQKwek5fG6tRYZkefudW4yvFf7jS+gE8pv0DEpN44AiX5OFdz59xTykbWNo53BfRhccR5StBUPuBuOvIowhAFX7ywhQ23sOCEh+hI8egD0DPaD9seGQgpFIJnvn1NH47cRPFZTJIlcDzvx+FoTwLqtDNUAYcRDTuQWTQ90hDAR7Lew9jZccQkbwJMBl8fA4vx68AkAVkXRqCEF9PNnailA0+xbCX2aqORelAizjgzHpWPLQabF6UTC4qHpJzy3A+nRUFPhPeQ37gC7hX1gp9YsSrH1oKiEqJeVtr0CKtJA0jOoaICogfxvyArLIsPLf/OZRYWDGqJLCd7ZiH1dpkDr4mo/uzsSG7XmMrXS7vU/NYCC7WxyccqCy3tXQ4yWvxr/Hb1m4KrqmWOxnWIQTDOlT9fa60sCi8+NsZ/PCg8EHypyM38PLGszZzObggyupUoqwQcZsE+5hrchSU6VGqreQFa1ykRtC0q6i8EiG14NXi4jeUcinkVVhs6hP3/xWJ8Mknn2DevHmYPXs2AGDFihXYsmULvvvuO7z44ov1vp7Zk39A2F8rUFSYjwnG3bx5nTHd0Y1KH5xuvwQ55YXQpL+L4yo5KiQSGAF4MUZ4GhmoGAb5MinOK5UoktkqUkbCYF36Y0AQUJ4yHZDqwKj9kGK8BlUIm78eVGlAd60Wf3t5YnJxCd7MyUMe441PKidil7E3psj2407FIYQwuZhQUoYB5TcxLJo1qys0p6DQnKrysy7IL8DDBeYv7tR55xDhF4KNXkqUaCvxyI8JOHAlB/9czMLK/dew7O9LKDX9x1y+/waAGKjkrTFz7iKEb5iAlensXSrdfzB+Mo5Bl9JD8DAUY0CUGv6jX0Sv8B5A+VOAWoMZEgm4bPq/nx6C2auO8gGLn97bHdpKI1YdTMIrt7PiAQDGdg3D7ydvIq9Uh4+mdrMJaJrRvyW6hPti0c5EpAI4oI4BwLY7niffivGyI9gcOBv/pHtgouwAhqadxqWP26Jzj4GQqr3ZAFGpHDnXLuMNGVtnYPSN4wDX4LP/fKDLFODYN8DptehnSMA3t/vjlp5dof5jHpC3jX27650Bk3hgfMIhuf8XYN19QNFNoP1Y9iZUng+snsjeNCyoUAqfuO6M64KXBvTjP+vMgS3x139pYAymm4QpiFAZcBAAcINZj3Pn2GDER/EUBhrOYapsH8oZJdpK09BXas4gCMnYx1tzeDpOEAZhxtwCZyiu0GPEJ3sBsOV3p/RjhVEfR+fozM2Q8srN/SjeiH8Dm69txrSO02zOifaJRs/QnjiYyn5ersukUwS0Eu7HDGbdFa5gXaMhNQGI6CU+1xm0pv9//jHAjI2ArHpFpLZd38Zvy6XCr/olg5fg8X8ex4t96//71FWUFmmelpkNf5/PhL+XkhcPgHhGh9mF4bqAsLZadIvUoG2IDxZP6ITFm8/za+oW5QfAtmplcS0FVJZpzRkYDYkGJyB0Oh0SEhKwaNEifkwqlWLkyJE4dEg8YK6u6RTeEZ0e+pTdMegBqZw18Ror2f/cEil6cybGY1L2ZpJlqggYezsQdzegiYS+MB3M8e9wLP0g8mRSFEml2OfpgUpIcEalRIWU/eP3iFojuo4P8ioQIG+LWYVnoFHE4J/h3yO4VXfM0qgxzcggxGee2f9Xmoug0+uw7Z/X8IWfHzb7mDscttXpsCIjG28EBeBfTw+MLinFC3n5MECCUJkHMOhJwDsU6DUbEUqzS8RbJceauf0w5MM9uJFbhne22vaXCPZRYcX0XmgT6Q88eQz9d7wEQAKMfgevSyQA5tsGFnn42VxH46nA748NQmpBOQrL9OgUzsZTWJq0AcBHrcDPDznuLdAj2h+fT5mAOzd9j0LPMlS+lIHiU3/Ad/uTCDfm4uG8j/CwRaXg8PKjQPxRwTVECxGPeR/o+xAglQKRvYCcS0DqcYz0TgJ+eRu4cdA8N8uUfjnxC0i63ceeM/8wcH4T0HkywP2eJ/4P+Otp801k1DvIkZcBF37kL9W9RSt4W5jo4yL9sP/5YVi4/2f8l5uKebcBWsN/2GB6UpdKpBjVKRR7E7PRr3UA2obcjqcPdjGdzWCy9ABGyhJQwnjiJNMWHWWpmNm2HGg/BghoXa0MjpwSLXq/be6wOahN1b5ahmEEFohifTG0Bi1UMhU2X2P7rFinbAKAziToOwZ2hEwiw/XC60guSka0b7TNXBtkCmDiF2xH0Lu+daphlQ2eVpaXlcOBF5JYN1R1qDD9DtS+1XJdAIDOopZEO39bK0vfFn1x6P5DNc7Wqg8sYxq4YMntZzPwyJoEm0JNYhYIPgujGjEQXGwFACy5syuGx7IWk1mDWuHX4zdxPr0Ix5Ly8PSvpzB/aFu+uBMH51qpKWUWNSAaEg1rNQBycnJgMBgQGir074SGhuLiRdte71qtFlqt2YdeVGS/G1utwPltVXbyA/rMYV8iKCIAdLodg4ozgMoKQFeKaWd/A/xbobIoFVtOfYXlPmqky23/WZ4tqkCfeXtZ/2ylDpArEeNonV6BwMAFiGw7Eu/Gf47XruzCZ0otDnmo8YaqNUJHPI4v/v0YyExmA9/63sd+4XW7l02/c8BjQ9vghd/M9QaUcin+evwW+Hko4OepND8xSCTAmCU257sSlRzh54EIP/GS1a7QWtMaXgovlOpLcST7BAb1vQ/oMBjY9wFwgs0WgWcQMiWB2FoYAwYSeKICWiiggAGRkmykMkGY2CUQnhc3AF2nAv2FxbMQ3oM1/f9hJ8B16Ets9D6Hyke4DwBd72Jfu98C8q4CfeYi+/AbgiktvFrYXDrUV42uIR3xX+4J/JS4SnBMKZPh6+m9UViuh69aDomEzRpZFZ8EQIKNxsGY8cgLePWPs2w3QwMw+e5RNepHEX81V7D/qEWgmT3ytfmC0uUAkJiXiJa+LR2e19KHPR6gDkD3kO5IyEzAzO0zsWnSJvgonbAfc1kV1UXEdYMru9l/x+rAiUeLWg2uUqAt4Le/HPGl6JzGIB4AQCEzf19wN/Qj19m/L+uW7GIWCA9F9V0YXOaHj1qO+/oKBam/F/v/4+0t7IPUM+tPIzZM+PdWUen6e4rBrd2dRaPEaHACwlWWLFmCN954o+qJDQmfMPN2KFt2WA5g4tAXMT75CDJP/4jyvGt4WncNEQYjPu40F15d7zEHd7lSFz8kFpi0HGqGwfN519jI+g7jALkKGFi9wkr39ImGj1qB7GItJnQLh7bSIBrY1JCQSWWY2GYi1l5ci/WX1mNQxCBWKN3xGdBzJts2ue0IhEikKN93FR/tSIR1/ZgvpvWEZ3sfIHEM0PF22zdpM1xYlbDrVDbKf8Mc9kn+lqcBsP0slp9ajjcHvokOAXYKBI0wdxm9UiB0+luXIuYYEzNG/AndoIPBaBA8rS2+ozMW39EZFzOKEBPoBbVChq8f6I1B7/0DAIhbvBNr5/XDwDauVSms0Bvw/cEkvL/dLPaPvjwCIT5VP0lzHUZDPUPBMAyyyrMwbes0PNrNccYR1/cEALoHswIipzwHP543Fz2rU7xFgtluxFdfQFhaIKpJfgXrpgtQByDUy/3BdjXB0qrA3dDt1YQQy3jj4hO0lWwmgzOp93qDEQ+uMsc4iVXL9POw/R4uMgkahUwCvYGBttYsEO4vWy1GgxMQQUFBkMlkyMwUlh3OzMxEWFiYzfxFixZh4cKF/H5RURGioqJs5jUW5NH9EBHdDwCwOf8G+yVSXVOoJRIJENiGfdUC47raPgU3dCa1nYS1F9did/JufH7yc8zvPp99Cos0+6slAB4b2haTe0TgUmYJ9l/Kxk9HbuCnuf3NhaPipoq/QZvhrA+9LA/oPZst4CSRAHN3CaYt2L0ADBg8uedJbJ9iG1Brid6gtxEQkT7iFqLuId3Rwb8DEvOFVREZMDiUfgi3RNjGLsSGmW9SEX4emD+sDZbvuQoA+GhHIn5/zHkBkZxbhpf/OIN/L5tjECZ1D3dKPACstQEAonyikFuRC5iMEV+eFn+CBtjyzAPCzS6se2PvxbdnvwUAnMmpp6qcmgjW7SdTse6eDbOBk2uAMe9VrwlWLVgg8rWsgPBX1cJ3h5uxjIHgik+V68RvzI5cGABb08HbieZT59KKBH/HYtf187S10KUVsvFHQd4qpBdWCFwgNYEsEE6iVCrRq1cv7N69G5MmTQIAGI1G7N69GwsW2D4xq1QqqFTOt7ptVPg7Nt0SrhEbEItwr3Cklabh6/++Rjv/dnaLbbXQeKCFxgND2gfjlfEdnXO7KNTArL+qnMaA/RJMLUnlx3Yk7YBapsaQqCGCuRmlGTY5+yoHrZ1X3LYC/zv5P0gkEhRqC5Fakorzuedxo+iGqICw5rnRsbiWXYptZzNwIb3YJl7FaGSwfM8VeKnkmD0ohj92KqUAk5abYz76tw6A3sDg6ducT2ncmcSWIB8UMQi7b+x26pxWfsIgyDCvMPw49kfM2DYDl/MvO/3eNeY2Uy2J0lw2LsqgZQtWTd/g+rVqwQKRWco+gLm7z0VtILX4++NKmFfYKe4kl9n+P1XJpZBI2NYtZdpKpwREmVXZa7H0UH9P++LQ31OJ9MIKu+t0FbMFomHdshvWakwsXLgQM2fORO/evdG3b18sW7YMpaWlfFYGQVQHiUSCDgEdkFbK9le4kHvBqWqdtVFJzhGX8y/j2X1soadlQ5dhRMsRANgvS26toZ6hGBMzBne0vcPudQD2hrF44GJ+/4NjH+B87nmn2qtzfHpvD+x8bTvK9QZkFWsRasp4YRgG89ee4CvtRfh7YHTnMBy6msvnxQPAzAEt8cbELqLXdkR6KZtK2zOkJw6n2RaIkklkkElk6BDQAc/3eR5rL6zFwt4LbeZxLp6c8hwYjAbIpPX41OYVyJb9/vcj4MouIPM8ENrJtWtwmSjKKuuw2mVvyl4AQOsaVF1tiBgZ9u/Q3o3ZKNLJVCKRIMBTidxSHbKKtXwGlyPyrSrDilkggrwdCAhTfERtB1G6s3W3GA1rNSbuueceZGdn47XXXkNGRga6d++O7du32wRWEoSrPNztYexJYds1cyW13UlqSSru3GTuwbDk6BKMaDkCeoMe9225j3dHtNa0xrN9nrV3GbtwHRjTStKqmGlGKZciOsAT13NKcTWrBKG+alzOLMaWM+mCMr0nkvMxunMYPtghDG7uLNJy2xm4wD8/lZ9oeeUJbSbgtQGvQS5hA0G7h3QXvU6AOgASSGBgDMjX5tf/U/iIV4HMs8Cl7cCed4B7beNSHMIVklJ6OZ5nh5zyHPydzGbAtPevYVGrBkBUgCcm94jAxpOsxU5babRbXtreE3qrIC/klupwNbtEtCW8NfllwgJRSpEYiBYOgrv9TNaJ2nZheNZCUarapMGG4S5YsAA3btyAVqvFkSNH0K9fP3cviWgCdA7sjB/GsFkX/2X/xwebcWSWZkJr0IqdKsq1gmtYc34NssuynZpvfe1VZ1cJ378sEylFKbhRdEMQyxDhI6xu6SwR3ux5XICis3DR5L+fTEWF3oAxn/6LZX8LXQLHk/JRUKbDSVOtDo5b2rp+w75RdINP4fRT+9lkYwCARqmBQqqo0iIkl8r5qq5ZZVkur6VW6PkA+/PyTsfzOIwGIOMs28BMb2oGVs0eFecsOrbe1rKKLqmNhI+ndgMX+1hUrrdrgbBXRbZ1MCvGLIs8OSLJap5YEKW97DCpBPBVs0Km9iwQpjoQKhIQBOFWeob2RIxvDABWRHAkFyVj5IaRmLdznlPXySnPwcQ/J+L9Y+/jf6f+59Q553PPC/Z/TvzZZs43Z7/hzfkcbTTVC35tpWFjBK4XXufLGjvDvaaUtQ0JNxH76nZBa/GFpriGhBv5fJnumEBPnHrtNux9dijCq5F2a9nbxVfpi7yKPJs59oJHxQj2YIs7ueK6qVVaDmJ/GnRsR9Kq2P0GsGIQsP8Dcy8MRfUymy4XsEJvbKux0KiqZw1qaEilEj52oaiiEuV2bsxRAeK/swAvNm6opKJS8Ldsj28set4A4hYISwEhKGujkPHZIOuOJlufVi1KteZOnA0JEhBEs6RbcDcAwOlsc0tmrnLfyayTfLCWI96IN6cPH0o7hGMZxxzMBtacX4MHtj0geuyX23/BsqHLAAC/X/4dj+0Wph/GaGKqXI8YkT6RUEgVqDBUILU4teoTTAxpH4xZA23fUymTYt7g1hjSnr1Bf7iDtZL0jgmAn6cSMUHVM7tbWkjkUrmogJjcbrLNmD3a+rF1J7jMjnrHsisnFxTpiIOmQnV7l7A1YoBqWyC4AMpIb+cFV2PA15SGXFShh1bEAjG4XRDmDhaP+eAsCLsuZKLz69uxIeGm6Lyb+WVY9vclm3ExAeHvpcTDQ1rDWyXH2C7mDEEPpZx3XaQXVtjUqqgOx2+w/x+8VdWvzVIXkIAgmiXdQlgBsfLMSpSZnvgs3Qu5Fbmi51lyvcj8lJJemo4HdzyIhMwEu/PfP/Y+v31Ph3sExzoFdsLAiIFQSIVfEHKpHE/2fBKDwgdVuR4x5FI5OgawlSTHbRyHlCLnXRmTewjdJhF+Hvh2Vm94KGXo11rYL2Jqr+rfrG4W236Z39vhXgBAj5AeGNdqHH4c+6PD7BNrOgWygYsbr2yE3lg75YRdQioDuCJWWicEhCVcDEQ1LRDc37F1C+/GDueeKK6otImB6BapwY9z+tnNsOAKTN3ILUOF3ohn15+2mZOYUYzbPz9g46oDxF0YALBobEf89/ooDLXo6eGhlCIpx9xRNTm3TOxUp6nQG3A2lf0bGtmx/ntBOYIEBNEs6Rtmbsiz7MQyfHHqC6w8Yy4CdTbnrNhpPFqDVvTGtyd5j2D/WMYx7EvZZzNvXKtxNmMecg908BcWlprYZiLmdp1bo0yQ/uH9+e0n9z7ptCsjtoWwqt7bk7tgcDvW8jCui7kOSISfR43aCk/YOIHf/nQY+yQ+t+tcrBi5AitGrsD7t75vN2DSHmNbjQXACrsr+VW076wrVNUUEJwFopoCosJ0viuCqzHgY4orEIuB8FY7zgcQKzBlScKNfIxeth8FZeJiU+ngfKlUIijS5qGQwcsiViEp17m4C3t8sstsEamuha+uIAFBNEta+rbkrQDrLq6zKVb0xakvbNwYq8+txvSt01GoLURSYRIMjAE+Ch9BqeWMMrPPvdJYiQd3PIgF/ywQNHhSy9ToHtIdz/Zmsyqe7/M8f+zuDncL3lOsbLWrPNjlQf69LudfFgTZOUIll+Hgi8MR6KXEmM5hGGwRHBkT5MWbdYfFBtu7RJXoDXpBa+kANWvZUMgUGBQxqNpP0YEegbybKrm4dvzQLsPVcSjOBJKPsAGSdrEQiJwFQl5NAWFgBYRYJktjxtfCAlFqVZbapwrTvrUFwcfKUrHtjDDmyBqxNE6xtQGsC+Pl8ebU3dQCkfbsLvD1/mv8ttyJKpr1CQkIotmysNdCG5cBx4W8C7hacFUw9uHxD3E6+zTWXljLB6q182+HTZM28U/Olj53y8yMNefNDdJ23LUDUokUD3R6AH9N/gvTO07nj01uNxnHp5vTS+2VrXYFL4UXZnaeiX5hbCaTpeulKiL8PHD8lZFYMaOXTRvh7U8OxoODWuG5UbHVXpt1dohTvSucJNqHDQTdf3M/76aqV7hKkuvuAb4bBZxYZX+uZcom78KongDgLBBqWVMTEOxNP7Oogi9pzRHgoCYDYCsgrHtmeFVRX6GqDtpCC4QUrYK8MHtQDADYtWpUh7quSeMqJCCIZounwhPzugozLj4c8iEfuT5502RcK2TVv+UN6N/Uf/HWobcAsHn2UokUXYO6AmDTEbkURO5cADiQegAA69vnnrIlEgla+ra0+VJQyVT8nH4tai99mbOUJBUmuXSevS+t1sHeeG1CJ2hESvo6i7V1oDYFRIgn6y/edHUTZu9wQxE66y6z8XYydc7+BujMbarBxeJU0wLBxUA0OQuE6SadkmcrBu+x6tJrjXUlSb3BbF1kGAaf7havWtqxBSsCB1WRmmyZkcFZIzhRURtBlA2VBllIiiDqCy7YDgAe7fYoxsSMwXP7nuPHjmccR2tNa9wousGPWfZYmNaR7eIY5BEEb4U3SvQlSC9Nh6fcE4/8be7WydV0cFYQbJy4EaW6Uv4mWBtwAsLys7gbyzoNPkof+Kn8au3alu4P6/TZesGyaR7ABlaKseFB8fFqWiA4AdvULBBcDESyiICIi3ScrmptgbBsxnU5q8R6Os/PD/VHuc6AMI3j36WliOZKb3MCoqgJCwiyQBDNmrjgOHiYnvT6hPUBYE7xBICbJWygpFhBoud6P8enV0okEt5yUaQtwpZrW0Tfz9nSwgHqAET51m5TOG6tlgLCldoQdQFXp2FAiwFYO24tlLJqNJ+yg4fVE/yP53+stWs7hY+V+ynnElCpE58rBlkgBHBP9pyA8FXLEeHngceHt63StG8tICotLBAlVn0vLNF4KKoUD9ZwTfc0FmmnTRUSEESzxl/tjz8n/ok/J/7JC4gPb/0QnnL26ZXLtLAMguQY11qYSeGrZM2dRboiZJZl2syXS+SI8nFfp1hLC4SRMeL9o+9j0LpBLrs0ahNOmPUJ61PtWhf24P4NOZYmLK3V61eJWBfdk6udOze4IyCrnoG4qcZAcFUYs4pZgRQb5ouDLw7HM6M6ODoNgG0dB71FQGttuRg2LRiEp0a2w0xT/ZTacGE4U4/GnZCAIJo9LbxbCCwDLbxb4P1b2ZoNu27swvBfh2N3sm13yEC1MHWRs0AUagtxMe+izfxo32i7QZv1QYR3BDzkHqgwVODPK39izYU1KNGXYP/N/W5bEycgQr1qv8+NdQZHbQSkukS0yV1laUm4YvV3ZLDz9Nv/EfFxJ2iqWRjWVRi5hlXOYJ3GaXlfLihzwSrkgLhIPzw1sj0vVmpDQDhTNdOdUAwEQYjA9ZAAgOzybGSn2va6sDabchaII+lHcDLrJAC2BPXVQjabw7rGQ30jl8rRLbgbDqcfxmvxr/HjORW21pX6grPU1GasB4e1CyOrLMumPXmdEt4DeOwI4BcNXN/PZmMUWVUDta4RMeodICcR6HZftd+Wt0A0MQHhoRDerkJ8nP98YoWguL+F/NK6cTF4mARPua56DbWuZpcgrYYpoHUNCQiCEMFSQFgS6hmKzLJMPN7jcZtjnAXiz6t/8mOdAjvxAsLVYkh1weCIwTicLmyVnV7iOAe+rohPjeczVepCQFi7MMory1FeWV6/FRpDTCmuXEBlsVVvjvwk4f7ABTV6uwOpB1BWWQapRAp/lYgLpRFjbYEI8HI+XkatsBUQBWV6+HspUWCyEEzpGYl+rQPw8c5EZBY531DPHpzVw17jL0eUaisxdcUh5JXWjnWkriAXBkGI4KnwRNegrjY3oWf7PIufxv2EuV3n2pxjnUHwaLdHBab5hiAgxrQaYzNm3XAqqTBJ0GSsrnj474f57VDP2ndhWFsgAKBUX7OqgNXG1+Q+Kck0d9sEgF2vic+vBslFyXxTsrigOHgrvWvt2g0BawFRXGE/+NEasUqUR5PY/hLcDT7QW4m7e0e5ZNlwBCdatJWuByr/ezmnwYsHgAQEQdjlm1HfYMeUHXjnlnf4sTDPMMQFx0Eqsf2vY/0U3cavDXqF9uL32/u3r7vFOkmIZ4jNjbVAWyDYn7drHqZtnSYa91FbWPen8FLUfoleldy2lHOJ3n7KXp3iGcS+AOD0OvO4rvbWcyr7FL/dzr9drV23oeBhJSCGdnC+AqqYC+PhHxNwMjmfL0rFVZt8bCjb+XZ8XM2qwKpNtSe0lUaXgyHzaykuo64hAUEQdvBUeMJP7YfWGnOAZZCH/YIy1gIi0icSg8IH4dnez+KzYZ9BLm0YHkNr07algCjVl/IWid036k5AFGoL6+zaHJHekQhQB6Clb0vewuE2C4RUCvQ2FbNKOWIet7RG1BBLYSgmcBs7nkrz/x8PhQyD2zku7mSJvV4YX+y9Cp2pJoTCJCDGdm2Bf58fhs/u7VGD1QpFi7NWiHKdASl5ZShxwbriTpreXxlB1DIxvjH8drCn/acef7X5xnxf7H3oFNAJEokEMzvPxLDoYXW5RJeYG8e6X7jqmYXaQhiMrBnXsvw2F81fFxRUFNTZtTk8FZ7Yduc2/Hr7r3yFS2csECnFKXh+//OimTTpJel4Yf8LOJfrXD8RAeGmG9LpdUDSQXa7xMJ9FHev69c0UaYvE6yXaybWlLB0YYzuHOpSMKx16WoOy7LYlqmeUQGekNWw74TaovqlVu+cgHjy55MY/MEe/HLc+a657qRhPBIRRAPGW+mNVWNWAXDc4bCDfwf4q/wR5hWGl/q9VE+rc50p7aYg2icasQGxuOXnW8CAQZGuCP5qf2SXmwVEXkVena0hX5vPb1u6iGobLmCSc5FYu2vEeHH/i/gv5z/sS9mHI9OOCI6N+m0UACA+LR7/3vsvDEYDZPYqTFoTZVGFdM+7wPTfgHLT7+GRA0BIZ+euI8LcnXP5CqlRPlEC11lTwdKFUVV3TWssm2FJJcCWJwZj7Kf/4kZuGVqZOlwqZLWbnaOQSSGTSmAwMqioNECDqtNOd55ns5KuiFTHtCyX3VAgCwRBOEGv0F5Vfil7Kjyxbco2/DT+p3paVfWQSqTo16IfNCoN/2SeX8HeyCwLYNWlgODcJHFBcbijzR119j4cqSVs+qRlmXJ7cGWvyyqFJZN3JO3gtwu0Bdh8dTP6r+3vfB0NryBg+u/s9o0DwCcd2W0PfyC0C+vmqCaW5dU7B1ZfiDRkvCxcGOUuZjZILawJRgaICWRFQ2G5HrklbLyBdbGp2oBzY1QnE8Oa3x4dWONr1DYkIAiiFvFSeLm1WJSrBHuwLpmscragk2WZ67oUEPFp8QBQb0/KYpVE7WEvVsW6FPZLB15ChaECbx9+2/mFtDCXSUe56fcb2gWoxdoUnFhqasikEjxqCnAc17VmAY4eShl/c88pYVM2q2rZXR0sAylriqsltesDEhAE0Yzhggvn7ZyHxfGLseL0Cv5YobbQJluitricz3Y/7B3Wu06ub83yEcv5ba5XhD0sBcSIX0dg4+WNAOwHYLrUQdQz0HYstGYWg0qjMOCulaZVja7XkHlhTCwSXhmJMV3Cqp5cBVxzLi7jQVEHAsJVC0RdWEHqksa1WoIgahXLzJHfLv9mc/yHcz/UyftyrpIwr5rfCJxhcMRgKKVs4SGxxmiWWDb0yirPwmvxr4FhGFwpuCI636US2WKWhhoKCEtXy5R2U/B0r6drdL2GTqC3/TgkV/BSmQSEqRJlXdy8OQtEhRNBlEYjwwd0NhZIQBBEM0asAuTMTjMRoA4AAHx64tMat8LOKc/Bwr0Lse36NgBsqWUumLG+BIREIuE/q2WmiTWl+lLRFNPV5+03wdqbsheX8i9Vf3HhPat/LoBSHWsZUUgVWDxwscNUY8KMt0lAWKdx1iauWCDE4joeGcK6bD67r2YppXUFCQiCaMZY38D7hvXFs32eFdxEXYkfEOORXY9g141deH7/88gtz0VaaRoAtm6Bj8IF838NCfBgRREXMCrG6ezTMDC2X+Rfnv6S3xaLkZiyaYrzC+loETQ6/FUgrIvz54rAuVbqohhXU4azQHCIFZuqKZybxJmqmWIC4plR7XHwxeG4o1s9N4JzEkrjJIhmjLUFgmv53b9FfxxMY2sVlFeaG/qU6cswc/tMRPtE4+OhH1d5/dzyXCTmJ/L7Lx14ia9R0N6/ff01toK51LijVM5LeawlIcY3BklFSfw4Fzfx7i3vIlAdCL1RDw+5B+bsnOP6Qu74HGgzHOg0EfAMcP18K7jaFiQgXMPbSkDUhQWC69eR50RlSbGmWwqZtEGmb3KQgCCIZox1S/JRMWydg0e7P8oLiCJdERiGQW5FLob9yhbEuph3EZXGSv5pPLM0EwmZCRjbaqxAFKSVpAmuH58Wz2dg9A6tnwBKDmsBcTzjOLsOUyBnia4EHyewoqhPWB+BgOACFUfFjOJrgVi7dgq1hXxDNYd4+JmrUtYCZIFwjagA9oZsKyBqX8zyAqLECQFRC6me9U29ujDeeecdDBw4EJ6envDz8xOdk5ycjPHjx8PT0xMhISF47rnnUFnZOMp6EkRjwzIAcPmI5ejfoj8AoFtwN0xsMxEAe2NcfX41Lx44Siz6ONy56U688O8L2HVjFwA2BTS3PJdPDxWjvosdcTf39ZfWY8LGCZi9YzZm75iNE5knAEAQJClWCtpL4SUoJBblEyU4fr3wel0sW5TPTnyGxfGLUWms5NNtubgVQhwfk2CYEMf+zVu7MOoiiNLfkxUQzvS2qG7bb3dSrxYInU6HqVOnYsCAAfj2229tjhsMBowfPx5hYWGIj49Heno6HnjgASgUCrz77rv1uVSCaBYEegTiu9HfQS1To2twV8Ex7oabXZaNtRfX2pxbrCuGn9oPKcUpKNIVAQB2J+/GsOhhmPTHJEgkEtGOmBw9Quo3MIyzQFjXSTiRdQI9Q3sK0js9FZ74+66/MXLDSH7MuoeIj9IHvUJ7ISEzAQCQVJRULx1XKyorsPLMSgBAXHAcH69CwZOO2Th/IP65mIVZA9k0V2+VsJplXbowcp3orHktx02N3mpAvVog3njjDTz99NPo2rWr6PGdO3fi/PnzWLNmDbp3746xY8firbfewvLly6HTNY7uZATR2OgT1sdGPACAr9IXAETFAwCM2zgOpfpSQfGprde34lDaIeRr85FXkcffrCe0niA496V+L9V7u2lrdw0Hl9bJiSCAzUQJ9QoV9EHhgjAt+WLEF7wr5mTWSZuaDHVBbkUuv/16/Ou4WnAVAAmIqmgb4oOHbm3DWxq8VcKCb3URRKnxYN+jqLzqeirHkuwH9zZUGlQWxqFDh9C1a1eEhobyY6NHj0ZRURHOnatG8xqCIKqNM/78cznnkFIsbPwzf/d8m3mzuszitz3kHrgv9r4ar89VrAssTW0/FQCw7uI6lOhK+Cf5IZFDEOjBio22fm35+QEqWwHhqfDEgPABAIDfL/+OR/9+tE7Wbkluea5g/8+rfwIgAeEqXlYWCMvmV7WFuQ5E1e6J3BLHBc4aIg0qiDIjI0MgHgDw+xkZGWKnQKvVQqs1/+KLiopE5xEE4RqcBcIRRbqiKn3/bf3aor1/e37fXW3NLduyA0Abvzb89qH0Q3jv6HsAhMLJMs5hXtw80etaBi8eTj9cK2t1xKcnPhUddyqAk+DhUiw5Ar2VdmZWH15AOFEgqqwRxkDU2ALx4osvQiKROHxdvGjbFre2WLJkCTQaDf+Kioqq+iSCIKpE7IYU7hWOEA9z6mehthBH048CYM3+YliKBwBu6xXip/YTiJfxrcbz26/Hv87HQFgKJ8unerGiW0D9Zj8wDIOjGUfdvo6mgGUQpVIuhaey9oWt2tRGXGtlgTAYGTy3/jTWHDa7/xpjEGWNBcQzzzyDCxcuOHy1bt266gsBCAsLQ2ZmpmCM2w8LE69Yt2jRIhQWFvKvlJTG0UedIBo6YhaI9Xes51ubA2xK5LXCawCAGZ1mYFDEIJtzJrebLNhv59+udhfqAm8OfJPf1qg0mNyWXVuxrpgft/zcarm5gZG9LAfrG/f2pO21slYx1l1cZ/eYp9yzzt63KWIpIOqqhLQ9F8buC5lYn3ATr/xxlh8rbYQCosaSKzg4GMHBwbWxFgwYMADvvPMOsrKyEBLCqv1du3bB19cXnTp1Ej1HpVJBpaqd2ugEQZjxVQkFxPuD34ev0he+Sl9M7zgday6sQXppOhgwANgn/DaaNjiYepA/p1NgJ/QL6wcAWDVmFVafW40X+r5Qfx/Cii5B5qqPluWtLbH83JZiwrJHhiXWAuK5fc9hQIsBdeJS2HR1k91jZIFwjbrovmmNWi7eC6PAIqhSbzBCIZOiXMcG4KoVUlTojbizR0Sdr6+m1KszMjk5GXl5eUhOTobBYMCpU6cAAG3btoW3tzdGjRqFTp06YcaMGfjggw+QkZGBV155BfPnzyeRQBD1jOXN85Fuj2Bc63E2x35J/IUfU8lUGBQxiO8bsffuvVDL1XxhqV6hveq99oM1rTSt8OPYH/kgSe6nJZafe3j0cAwKH4Ruwd1s5nGI3bgv5F3ga2rUFmX6MlzIuwAAeDjuYXz131eC445SZglbgn3M95T5w9o4mFl9OBdGRaXQumA0Mvx2bokOYRo1HwPx/ay+yCquwMiOwnjAhki9CojXXnsNP/xg7u7XowebB75nzx4MHToUMpkMf/31Fx599FEMGDAAXl5emDlzJt588017lyQIoo7wV/tj8YDFUMqUmNBGmIYZ7CludRwYPhAfDfkI7fzaid6cGwKWtRo0SlsrgaXlQClTYsVtK2zmWOIltxUQiXmJtS4gMsoyYGSM8FZ4i7qByALhGu1DffC/+3ughcYDvVr6V31CNbDnwsgvM1sgsou1AgERplFjQJuG+X/HmnoVEKtWrcKqVasczmnZsiW2bt1aPwsiCMIhU9qLN4myzmiwZHTM6LpaTq0j5mZw9UYsVs/Csgx2Tfkn+R+czTnLF48K8Qzh3UaWeCooBsJVbo+r2yZVKs4CoTeCYRjeGmeZsplTogXDMCgzuTA8lbWfTlpXNKg0ToIgGgeWKZCNGTEBIVbG2hGhnqGY3nE6POQeiNHE4OUDLwuKa9WUJ/c8KdgP9ggWLYpFFoiGh2VtCW2lkd8/mVLAj5doK1GsrQTn1fBoRAKiQRWSIgiicaBRafBw3MPuXkaNEXNhtPNzLUtEIpHghb4v4ImeT6CVL1us6ljGMUEX0+piMNpG5nsrvdE7tDceinsIj3V7jB9XSmu/jgFRM7ggSgDQmgIpy3SVSLhhrjpZrjdg82m26VywjwredZBOWlc0npUSBNGgWNBjAQaED8CHxz60W2SpoWOZceGl8MK2O7fVqMR2bGAsfJW+KNIV4Yl/nsDKUStrtD4xEfJQ3EOQSCR4vMfjAIAW3i3gq/St19bohHMoZBJIJYCRYQMpNVDgRm6ZYE6F3oDrOWxH1Tu6hUMqbTz/jmSBIAii2vQK7YWfb/8ZI6JHuHsp1cIy4yLSOxL+6poF0ymkCszpOgcAW5UypahmdWmsBcR7g99Dp0BhSvuktpMwPHp4jd6HqBskEolNIOWN3FLBnHKdAZlFFQCAKP/GlUlDAoIgiGaLRCLBq/1fhUKqqLX6FA92eZBvsPX7ld9rdK2ySuHTqmVvDqJxYBYQrAvjRHKB4Hi53oD0QlZAhGnUaEyQgCAIollzd4e7cWLGCfQJ61Nr17w18lYAwG+XfqvRdcr0QgFh3RCMaPio5VwmhgFnUwvx9X62cqtCxroqyvUGZBWxWRkhviQgCIIgmjVc3YwCbQEYxjbl0lksLRAv93vZbjVMouFi6cJYfSiJHx8RyxaKKtcZUGpK4fRVN66wRBIQBEEQtQzXl4IBgwpDhcvnr72wFgt2L0B+BRut3zGgI+6NvbdW10jUDyqLjpzZxaylYc4trRAXxWYAlesMfBGpumgpXpc0LrlDEATRCLBswlWmL3O5zPSSo0sAAHkVeQCoTHVjhi9nrTcLhe5RfsgxFZMq1VXyzbzqoiNoXUIWCIIgiFpGKpHyN/2yyjIUagurdZ0zOWcAiHdGJRoH5oZaBpSbMjG8VDJ4mKwNeaU6fm5jqkIJkIAgCIKoEzgBseTIEtzy8y3Yf3O/U+eJxUx0Dupcq2sj6g/OAqHVG3kLhIdCzlec5NwaEgmgkjeuW3LjWi1BEEQjgRMQ/6b+CwB4/+j7Tp0nVjwqLiiu9hZG1Ct8EGWlAWVac78LX7UCAJBlEhAeClmjKwbWuBwuBEEQjQTruAVne2yU6EtsxmI0MbWxJMINWGZhlJlcGJ5KGfQG9vZbXNH4mmhxkIAgCIKoA7hMDFcp0dkKiDCvsJouh3ATaouOnJwLw1Ml55tncTSmJlocJCAIgiDqAA9F9TInOAuEh9wD/ip/9Ajt4XKHUKLhwFkgSrQW2RYKGaydFZ6Kxnc7bnwrJgiCaAT4q4R9NZz1bxfrigEA0T7RWD9hfaPzixNCfEyxDlwFSoC1Nshlwn9XNVkgCIIgCAAIUAcI9iU2z5ziFGgLAAB+Kj8SD00AseqSKrkUSpnQquTZyIpIAZSFQRAEUSdYC4hrhddgZIxVnscLCLVfHayKqG80HgrB/hPD20Iikdi07W6MQZQkIAiCIOoAT4VtEOXelL1VnmdpgSAaP74WAiJco8aTI9vz+48ObcNvN8YgShIQBEEQdUBL35Y2Y8cyjlV5Htf/QqPS1PqaiPqHq/cAAH1bBUBmYXmI9DcH2nqQC4MgCIIAgFsibsEr/V7Bn5P+xBM9ngAA3Cy+WeV5XNlrskA0DXw9zDEQ1u4Mb5X5WGN0YVAQJUEQRB0glUhxT+w9AIBLvpcAAMX64irP49I4fZQ+dbc4ot4I9lHx2x5WzbIsrRPWxxoDZIEgCIKoY3wVbDMsLkXTmozSDHx56kvklOegTF8GAPBSeNXb+oi6I9jbLCAKy3WCY95qskAQBEEQDvBWegMQrzIJAM/uexans0/jeOZxlOpLAVS/kiXRsLBMxbVO3fRp5AKCLBAEQRB1DOeOsGeBOJ19GgBwNOMoyirJAtHU+HhqN/SM9sNjw9oKxgO8lPy23mDbhbWhQwKCIAiijuEERIm+pMpaEJwFwroZF9F4mdIrEr8/NgihvmrBeIiPGiM7hgAA+rbyFzu1QVNvAiIpKQlz5sxBq1at4OHhgTZt2uD111+HTif0Cf33338YPHgw1Go1oqKi8MEHH9TXEgmCIOoETkAwYJBUmGRzXC4xm7IpBqJ58dWM3oh/cTh6tQyoenIDo94ExMWLF2E0GvHVV1/h3LlzWLp0KVasWIGXXnqJn1NUVIRRo0ahZcuWSEhIwIcffojFixfj66+/rq9lEgRB1DoqmQo9Q3oCAHbd2CU4tvnqZlQylfx+haECAAmI5oJMKkG4X+O0NtVbEOWYMWMwZswYfr9169ZITEzEl19+iY8++ggA8NNPP0Gn0+G7776DUqlE586dcerUKXzyySd46KGH6mupBEEQtU6fsD44kXUC2eXZgvGXDrwkOl+skiVBNCTcGgNRWFiIgACz2ebQoUO49dZboVSaA0tGjx6NxMRE5Ofni15Dq9WiqKhI8CIIgmho+KtZH3deRV6Vc0M8Q6CUKqucRxDuxG0C4sqVK/j888/x8MMP82MZGRkIDQ0VzOP2MzIyRK+zZMkSaDQa/hUVFVV3iyYIgqgmXHMtrlQ1AOiNetG5z/Z+ljpxEg2eGguIF198ERKJxOHr4sWLgnNSU1MxZswYTJ06FfPmzavR+y9atAiFhYX8KyUlpUbXIwiCqAs4C8TxzOMoqCgAAOSU5Qjm7JiyAz+N+wljW42t7+URhMvUOAbimWeewaxZsxzOad26Nb+dlpaGYcOGYeDAgTbBkWFhYcjMzBSMcfthYWGi11apVFCpVKLHCIIgGgqR3pH89jdnvsGzfZ7F8lPLBXPCvcMR7h1e30sjiGpRYwERHByM4OBgp+ampqZi2LBh6NWrF77//ntIpUIDyIABA/Dyyy9Dr9dDoWBrhO/atQsdOnSAv3/jy5ElCILgiPSJxKS2k/DHlT9wMvskACCrLMvNqyKI6lNvMRCpqakYOnQooqOj8dFHHyE7OxsZGRmC2Ib7778fSqUSc+bMwblz5/DLL7/g008/xcKFC+trmQRBEHXGlHZTAADZZdlgGAZppWluXhFBVJ96S+PctWsXrly5gitXriAyMlJwjGHYEp4ajQY7d+7E/Pnz0atXLwQFBeG1116jFE6CIJoEwZ6stTa9NB2P/v0o0kvSAQBxQXF4sueT7lwaQbiMhOHu3k2EoqIiaDQaFBYWwtfX193LIQiC4NEatOi9prdgTC6V4+i0o1BIFXbOIoj6w5V7KPXCIAiCqCdUMtuA73Z+7Ug8EI0SEhAEQRBupFNgJ3cvgSCqBQkIgiAIN9IxoKO7l0AQ1YIEBEEQhBvpEtzF3UsgiGpBAoIgCKIe6R1qDqK8p8M96BRALgyicVJvaZwEQRAE8OGQD7E+cT3ubHcnQr1Cqz6BIBooJCAIgiDqkSCPIDza/VF3L4Mgagy5MAiCIAiCcBkSEARBEARBuAwJCIIgCIIgXIYEBEEQBEEQLkMCgiAIgiAIlyEBQRAEQRCEy5CAIAiCIAjCZUhAEARBEAThMiQgCIIgCIJwmSZXiZJhGABAUVGRm1dCEARBEI0L7t7J3Usd0eQERHFxMQAgKirKzSshCIIgiMZJcXExNBqNwzkSxhmZ0YgwGo1IS0uDj48PJBJJrVyzqKgIUVFRSElJga+vb61c093QZ2oc0GdqHDS1z9TUPg9An8lZGIZBcXExwsPDIZU6jnJochYIqVSKyMjIOrm2r69vk/nD46DP1Digz9Q4aGqfqal9HoA+kzNUZXngoCBKgiAIgiBchgQEQRAEQRAuQwLCCVQqFV5//XWoVCp3L6XWoM/UOKDP1Dhoap+pqX0egD5TXdDkgigJgiAIgqh7yAJBEARBEITLkIAgCIIgCMJlSEAQBEEQBOEyJCAIgiAIgnCZZisgli9fjpiYGKjVavTr1w9Hjx51OH/9+vWIjY2FWq1G165dsXXrVsFxhmHw2muvoUWLFvDw8MDIkSNx+fLluvwINrjymVauXInBgwfD398f/v7++H975x0eRbm28Xv7pvcKIaH3jkSQKqGrYANsgEdAUT7FjuccRSwHVBQVUWwIKqKoCCpKJ4CAIL2HFkJNQkJ6z+58f8xO3dnN7mY3u0me33Xl2pl33p19J7M7c8/zPiUlJcWq/+TJk6FSqSR/I0aM8PRh8DhzPEuXLrUaq9FolPSpb+do0KBBVsekUqkwevRovo+3z9H27dtx++23Iz4+HiqVCqtXr67xPampqejRowcMBgNatWqFpUuXWvVx9vfpTpw9plWrVmHo0KGIiopCcHAw+vTpg/Xr10v6vPrqq1bnqV27dh48CinOHlNqaqridy8zM1PSrz6dJ6XfikqlQseOHfk+3jxPc+fOxU033YSgoCBER0dj7NixSEtLq/F93rw3NUoB8cMPP+CZZ57B7NmzceDAAXTt2hXDhw9Hdna2Yv9du3bhvvvuwyOPPIKDBw9i7NixGDt2LI4dO8b3efvtt/Hhhx9i8eLF2LNnDwICAjB8+HCUl5f75DGlpqbivvvuw9atW7F7924kJCRg2LBhuHLliqTfiBEjcO3aNf5vxYoVdXE4Th8PwGZjE481IyNDsr2+naNVq1ZJjufYsWPQaDS49957Jf28dY4AoKSkBF27dsWiRYsc6p+eno7Ro0dj8ODBOHToEGbOnIkpU6ZIbriunHt34uwxbd++HUOHDsUff/yB/fv3Y/Dgwbj99ttx8OBBSb+OHTtKztNff/3lieEr4uwxcaSlpUnGHB0dzW+rb+fpgw8+kBzLpUuXEB4ebvV78tZ52rZtG5544gn8/fff2LhxI6qqqjBs2DCUlJTYfI/X701MI6R3797ME088wa+bTCYmPj6emTt3rmL/cePGMaNHj5a0JScnM48++ijDMAxjNpuZ2NhY5p133uG35+fnMwaDgVmxYoUHjsAaZ49JTnV1NRMUFMQsW7aMb5s0aRIzZswYdw/VIZw9nq+++ooJCQmxub+GcI4WLFjABAUFMcXFxXybN8+RHADML7/8YrfPCy+8wHTs2FHSNn78eGb48OH8em3/T+7EkWNSokOHDsycOXP49dmzZzNdu3Z138BqgSPHtHXrVgYAk5eXZ7NPfT9Pv/zyC6NSqZgLFy7wbb50nrKzsxkAzLZt22z28fa9qdFZICorK7F//36kpKTwbWq1GikpKdi9e7fie3bv3i3pDwDDhw/n+6enpyMzM1PSJyQkBMnJyTb36U5cOSY5paWlqKqqQnh4uKQ9NTUV0dHRaNu2LaZPn47c3Fy3jl0JV4+nuLgYiYmJSEhIwJgxY3D8+HF+W0M4R19++SUmTJiAgIAASbs3zpGr1PRbcsf/yduYzWYUFRVZ/ZbOnDmD+Ph4tGjRAg888AAuXrzopRE6Trdu3RAXF4ehQ4di586dfHtDOE9ffvklUlJSkJiYKGn3lfNUUFAAAFbfIzHevjc1OgGRk5MDk8mEmJgYSXtMTIzV/B5HZmam3f7cqzP7dCeuHJOcF198EfHx8ZIv2ogRI/D1119j8+bNeOutt7Bt2zaMHDkSJpPJreOX48rxtG3bFkuWLMGaNWvw7bffwmw2o2/fvrh8+TKA+n+O9u7di2PHjmHKlCmSdm+dI1ex9VsqLCxEWVmZW77L3mb+/PkoLi7GuHHj+Lbk5GQsXboU69atwyeffIL09HT0798fRUVFXhypbeLi4rB48WL8/PPP+Pnnn5GQkIBBgwbhwIEDANxzzfEmV69exZ9//mn1e/KV82Q2mzFz5kzccsst6NSpk81+3r43NbhqnITzzJs3D99//z1SU1MljocTJkzglzt37owuXbqgZcuWSE1NxZAhQ7wxVJv06dMHffr04df79u2L9u3b49NPP8Xrr7/uxZG5hy+//BKdO3dG7969Je316Rw1Br777jvMmTMHa9askfgLjBw5kl/u0qULkpOTkZiYiJUrV+KRRx7xxlDt0rZtW7Rt25Zf79u3L86dO4cFCxbgm2++8eLI3MOyZcsQGhqKsWPHStp95Tw98cQTOHbsWJ36ybhCo7NAREZGQqPRICsrS9KelZWF2NhYxffExsba7c+9OrNPd+LKMXHMnz8f8+bNw4YNG9ClSxe7fVu0aIHIyEicPXu21mO2R22Oh0On06F79+78WOvzOSopKcH333/v0AWsrs6Rq9j6LQUHB8PPz88t595bfP/995gyZQpWrlxpZVaWExoaijZt2vjseVKid+/e/Hjr83liGAZLlizBQw89BL1eb7evN87TjBkz8Pvvv2Pr1q1o2rSp3b7evjc1OgGh1+vRs2dPbN68mW8zm83YvHmz5AlWTJ8+fST9AWDjxo18/+bNmyM2NlbSp7CwEHv27LG5T3fiyjEBrHfu66+/jnXr1qFXr141fs7ly5eRm5uLuLg4t4zbFq4ejxiTyYSjR4/yY62v5whgw7QqKirw4IMP1vg5dXWOXKWm35I7zr03WLFiBR5++GGsWLFCEmZri+LiYpw7d85nz5MShw4d4sdbX88TwEY7nD171iFBXpfniWEYzJgxA7/88gu2bNmC5s2b1/ger9+bau2GWQ/5/vvvGYPBwCxdupQ5ceIEM23aNCY0NJTJzMxkGIZhHnroIWbWrFl8/507dzJarZaZP38+c/LkSWb27NmMTqdjjh49yveZN28eExoayqxZs4Y5cuQIM2bMGKZ58+ZMWVmZTx7TvHnzGL1ez/z000/MtWvX+L+ioiKGYRimqKiIee6555jdu3cz6enpzKZNm5gePXowrVu3ZsrLy33ueObMmcOsX7+eOXfuHLN//35mwoQJjNFoZI4fPy455vp0jjj69evHjB8/3qrd2+eIG8PBgweZgwcPMgCY9957jzl48CCTkZHBMAzDzJo1i3nooYf4/ufPn2f8/f2Z559/njl58iSzaNEiRqPRMOvWreP71PR/8rVjWr58OaPVaplFixZJfkv5+fl8n2effZZJTU1l0tPTmZ07dzIpKSlMZGQkk52d7ZPHtGDBAmb16tXMmTNnmKNHjzJPPfUUo1armU2bNvF96tt54njwwQeZ5ORkxX168zxNnz6dCQkJYVJTUyXfo9LSUr6Pr92bGqWAYBiGWbhwIdOsWTNGr9czvXv3Zv7++29+28CBA5lJkyZJ+q9cuZJp06YNo9frmY4dOzJr166VbDebzczLL7/MxMTEMAaDgRkyZAiTlpZWF4fC48wxJSYmMgCs/mbPns0wDMOUlpYyw4YNY6KiohidTsckJiYyU6dOrbOLg7PHM3PmTL5vTEwMM2rUKObAgQOS/dW3c8QwDHPq1CkGALNhwwarffnCOeLC/eR/3HFMmjSJGThwoNV7unXrxuj1eqZFixbMV199ZbVfe/8nT+PsMQ0cONBuf4ZhQ1Xj4uIYvV7PNGnShBk/fjxz9uxZnz2mt956i2nZsiVjNBqZ8PBwZtCgQcyWLVus9lufzhPDsCGMfn5+zGeffaa4T2+eJ6VjASD5ffjavYnKeRMEQRAE4TSNzgeCIAiCIIjaQwKCIAiCIAinIQFBEARBEITTkIAgCIIgCMJpSEAQBEEQBOE0JCAIgiAIgnAaEhAEQRAEQTgNCQiCIAiCIJyGBARBEARBEE5DAoIgCIIgCKchAUEQBEEQhNOQgCAIgiAIwmlIQBAE4TCpqalQqVRITU319lC8QlJSEiZPnuztYRCET0ACgiA8yNKlS6FSqSR/0dHRGDx4MP78809vD49QYNeuXXj11VeRn5/v7aEQhE+j9fYACKIx8Nprr6F58+ZgGAZZWVlYunQpRo0ahd9++w233Xabt4fnMAMGDEBZWRn0er23h+Ixdu3ahTlz5mDy5MkIDQ2VbEtLS4NaTc9dBAGQgCCIOmHkyJHo1asXv/7II48gJiYGK1assCsgqqurYTabfeaGrVarYTQavT0Mr2EwGLw9BILwGUhKE4QXCA0NhZ+fH7RaQcNfuHABKpUK8+fPx/vvv4+WLVvCYDDgxIkTAIAtW7agf//+CAgIQGhoKMaMGYOTJ0/y7z9y5AhUKhV+/fVXvm3//v1QqVTo0aOH5PNHjhyJ5ORkfj0pKQm33XYb/vrrL/Tu3RtGoxEtWrTA119/LXmfkg/Ejh07cO+996JZs2YwGAxISEjA008/jbKyMsl7J0+ejMDAQFy5cgVjx45FYGAgoqKi8Nxzz8FkMtX4P+PGmJqail69esHPzw+dO3fmx7Jq1Sp07twZRqMRPXv2xMGDB632UdP/8NVXX8Xzzz8PAGjevDk/7XThwgV+DHIfiPPnz+Pee+9FeHg4/P39cfPNN2Pt2rWK/7eVK1fizTffRNOmTWE0GjFkyBCcPXu2xmMnCF+ELBAEUQcUFBQgJycHDMMgOzsbCxcuRHFxMR588EGrvl999RXKy8sxbdo0GAwGhIeHY9OmTRg5ciRatGiBV199FWVlZVi4cCFuueUWHDhwAElJSejUqRNCQ0Oxfft23HHHHQDYm7tarcbhw4dRWFiI4OBgmM1m7Nq1C9OmTZN87tmzZ3HPPffgkUcewaRJk7BkyRJMnjwZPXv2RMeOHW0e248//ojS0lJMnz4dERER2Lt3LxYuXIjLly/jxx9/lPQ1mUwYPnw4kpOTMX/+fGzatAnvvvsuWrZsienTp9f4fzx79izuv/9+PProo3jwwQcxf/583H777Vi8eDH+/e9/4/HHHwcAzJ07F+PGjZNMOTjyP7zrrrtw+vRprFixAgsWLEBkZCQAICoqSnE8WVlZ6Nu3L0pLS/Hkk08iIiICy5Ytwx133IGffvoJd955p6T/vHnzoFar8dxzz6GgoABvv/02HnjgAezZs6fGYycIn4MhCMJjfPXVVwwAqz+DwcAsXbpU0jc9PZ0BwAQHBzPZ2dmSbd26dWOio6OZ3Nxcvu3w4cOMWq1mJk6cyLeNHj2a6d27N79+1113MXfddRej0WiYP//8k2EYhjlw4AADgFmzZg3fLzExkQHAbN++nW/Lzs5mDAYD8+yzz/JtW7duZQAwW7du5dtKS0utjnvu3LmMSqViMjIy+LZJkyYxAJjXXntN0rd79+5Mz549lf+BIrgx7tq1i29bv349A4Dx8/OTfNann35qNU5H/4fvvPMOA4BJT09XHMOkSZP49ZkzZzIAmB07dvBtRUVFTPPmzZmkpCTGZDIxDCP839q3b89UVFTwfT/44AMGAHP06NEaj58gfA2awiCIOmDRokXYuHEjNm7ciG+//RaDBw/GlClTsGrVKqu+d999t+SJ99q1azh06BAmT56M8PBwvr1Lly4YOnQo/vjjD76tf//+OHDgAEpKSgAAf/31F0aNGoVu3bphx44dAFirhEqlQr9+/SSf26FDB/Tv359fj4qKQtu2bXH+/Hm7x+bn58cvl5SUICcnB3379gXDMIrTCI899phkvX///jV+hniMffr04de5aZhbb70VzZo1s2rn9uvM/9AZ/vjjD/Tu3VvyvwwMDMS0adNw4cIFfvqJ4+GHH5b4s3D/b0ePnyB8CRIQBFEH9O7dGykpKUhJScEDDzyAtWvXokOHDpgxYwYqKyslfZs3by5Zz8jIAAC0bdvWar/t27dHTk4OLxj69++P6upq7N69G2lpacjOzkb//v0xYMAAiYDo0KGD5EYKQHID5ggLC0NeXp7dY7t48SJ/Y+b8GgYOHAiAnboRYzQaraYDHPkMW2MMCQkBACQkJCi2c/t15n/oDBkZGTb3Kf5cW+MPCwuTjJMg6hMkIAjCC6jVagwePBjXrl3DmTNnJNvET/TO0qtXLxiNRmzfvh07duxAdHQ02rRpg/79+2Pv3r2oqKjAjh07JJYGDo1Go7hPhmFsfp7JZMLQoUOxdu1avPjii1i9ejU2btyIpUuXAgDMZrNDn+Eott7vyti9QX0ZJ0E4AjlREoSXqK6uBgAUFxfb7ZeYmAiAzUEg59SpU4iMjERAQAAAQK/Xo3fv3tixYweaNWvGC4X+/fujoqICy5cvR1ZWFgYMGOCWYzh69ChOnz6NZcuWYeLEiXz7xo0b3bJ/d+HM/1ClUjm1X1v7FH8uQTREyAJBEF6gqqoKGzZsgF6v583dtoiLi0O3bt2wbNkySXbEY8eOYcOGDRg1apSkf//+/bFnzx5s3bqVFxCRkZFo37493nrrLb6PO+CeqMVP0AzD4IMPPnDL/t2FM/9DTkg4koly1KhR2Lt3L3bv3s23lZSU4LPPPkNSUhI6dOjgtmMgCF+DLBAEUQf8+eef/FNpdnY2vvvuO5w5cwazZs1CcHBwje9/5513MHLkSPTp0wePPPIIH4IYEhKCV199VdK3f//+ePPNN3Hp0iWJUBgwYAA+/fRTJCUloWnTpm45rnbt2qFly5Z47rnncOXKFQQHB+Pnn3/2yTl9R/+HPXv2BAD85z//wYQJE6DT6XD77bfzwkLMrFmzsGLFCowcORJPPvkkwsPDsWzZMqSnp+Pnn3+mrJVEg4YEBEHUAa+88gq/bDQa0a5dO3zyySd49NFHHXp/SkoK1q1bh9mzZ+OVV16BTqfDwIED8dZbb1k5Xfbt2xcajQb+/v7o2rUr396/f398+umnbrM+AIBOp8Nvv/2GJ598EnPnzoXRaMSdd96JGTNmSD7bF3D0f3jTTTfh9ddfx+LFi7Fu3TqYzWakp6crCoiYmBjs2rULL774IhYuXIjy8nJ06dIFv/32G0aPHl2Xh0cQdY6KIe8dgiAIgiCchOxrBEEQBEE4DQkIgiAIgiCchgQEQRAEQRBOQwKCIAiCIAinIQFBEARBEITTkIAgCIIgCMJpGlweCLPZjKtXryIoKMiplLQEQRAE0dhhGAZFRUWIj4+vMRFagxMQV69etarMRxAEQRCE41y6dKnGjLUNTkAEBQUBYA/ekRTBBEEQBEGwFBYWIiEhgb+X2qPBCQhu2iI4OJgEBEEQBEG4gCMuAORESRAEQRCE05CAIAiCIAjCaUhAEARBEAThNCQgCIIgCIJwGhIQBEEQBEE4DQkIgiAIgiCchgQEQRAEQRBOQwKCIAiCIAinIQFBEARBEL4Ow3h7BFaQgCAIgiAIX+bgcmBeMyB9h7dHIoEEBEEQBEH4MmseByoKgWW3AWazt0fDQwKCIAiCIOoL+5d4ewQ8JCAIgiAIwlcxVUvX0/70zjgUIAFBEARBEL7K0ZXS9Yoi74xDARIQBEEQBOGr5JyRrpfmemccCpCAIAiCIAhfpTibfe08jn0tyfHeWGSQgCAIgiAIX6XEIiCi27Gv5fnWVgkvQQKCIAiCIHwVzgIR2VZoW/OEd8Yiw6MCYvv27bj99tsRHx8PlUqF1atX1/ie1NRU9OjRAwaDAa1atcLSpUs9OUSCIAiC8F04AREcD4Q0Y5cv7fHeeER4VECUlJSga9euWLRokUP909PTMXr0aAwePBiHDh3CzJkzMWXKFKxfv96TwyQIgiAI38NsBkqus8uB0cDUzdJtXkbryZ2PHDkSI0eOdLj/4sWL0bx5c7z77rsAgPbt2+Ovv/7CggULMHz4cE8NkyAIgiB8B4Zh//Z8Apir2LaAKMBsEvpUFgPGYO+Mz4JHBYSz7N69GykpKZK24cOHY+bMmTbfU1FRgYqKCn69sLDQU8MjCIIgCM+z6VVg31dARYHQpjWwokKlARiTTwgIn3KizMzMRExMjKQtJiYGhYWFKCsrU3zP3LlzERISwv8lJCTUxVAJgiAIwv2U5AA735eKBw6VCjAEsssVxcCfs4CVk4DMY3U6RA6fEhCu8NJLL6GgoID/u3TpkreHRBAEQRCukXvWum3gLGFZH8S+VhYB5zYDJ1YDZXl1MjQ5PjWFERsbi6ysLElbVlYWgoOD4efnp/geg8EAg8FQF8MjCIIgCM9SnGXd1m+msGywCIiKIqCqnF3WKd8fPY1PWSD69OmDzZs3S9o2btyIPn36eGlEBEEQBFGHcGGbYsQCQTyFUW2Z2tcaPT8uBTwqIIqLi3Ho0CEcOnQIABumeejQIVy8eBEAO/0wceJEvv9jjz2G8+fP44UXXsCpU6fw8ccfY+XKlXj66ac9OUyCIAiC8A2KMu1v11sERGVxw7ZA7Nu3D927d0f37t0BAM888wy6d++OV155BQBw7do1XkwAQPPmzbF27Vps3LgRXbt2xbvvvosvvviCQjgJgiCIxsH1U/a3i6cwvGyB8KgPxKBBg8AwjM3tSlkmBw0ahIMHD3pwVARBEPULhmGgUqm8PQzC0zAMkLFL2hbcRLrOCYiyfMBczS43RAsEQRAEUTu2nMpCrzc2YcspBec6ouFw5EfgwDKg7Ia0/cGfpevcFMbWN4W2hugDQRAEQdSOfy3dh9ySSvxr6T5vD4XwFGV5wKopwG9PWW+Lbi9d5ywQEFn3SUAQBEEQRCPEmTwOXBSGGLV3buUkIAiCIAjCm5TlO95XryAgvIRPJZIiCIIgBK7mCyn8tWoVqk1mmBlAr6VnvwZBZQlwZgNwTObn0GEMENEaaDvK+j0G79a/EEMCgiAIwkf44+g1nMkqxpNDWuGbvzPwyprjku3DFmwHA+Cnx/ogIpAy8NZ71v8b2L/Uuj0wFhjysvJ7AiI8OiRnIAFBEAThIzy+/AAAILlFOH45eEWyrdrM4HxOCQDgj2OZeOjmxDofH+FGLvylLB4AwC/U9vsS+0nXWw9z14ichgQEQRCEj5FVWI6KKrPN7cXl1XU4GsIjLB1te5sx1PY2nRF4aDWwexHQczIJCIIgiMaOySyE5ZkZBldE/g9yKqpNdTGkGimrNEGlAow6jbeH0rCwZ4EAgJaD2T8vQ544BEEQPkBppWBVKCqvRkFZlc2+FdW2rRPuoLLajGlf78MXO87b7PPdnoto/8o63LN4l92Mw4QCNUVdaPR1MozaQgKCIAjCByirFKwKF3JK7fat9LCA2HQyCxtOZOGNtSclkSBi/v3LUQDAsSuFuJBrf7yEjNxz1m2BsUCXCWyUReItdT8mFyABQRAE4QOUigREWlYhAEClAvwUpgc8PYVRLZpO6TtvC3KLKyTbxdMtAPDt3xlu++zSymo8+MUeLNp61m379Dku7rZuM4YAYz8BXkgHguPqfkwuQAKCIAjCBygRTWGkZRYDAIa0i8HRV4dhaIcYSV97DpbuQKuWFu7alyHNlCgeKwBsOZXtts/+aucF/HU2B++sT3PbPn2OK/ut2zrfw2aU1NQf10QSEARBED6AeAojx/LE3zTMD1qNGvEh0loHnvaBKK6QCoQcmQVCHgVyJb8MZrN7/CD2ZziR1rm+kn1Sut71PuCWmV4ZSm2oP1KHIAiiASOewuBoEsqWaY4PlZZr9rQPhFwgXLoh+EHsu3ADR68U8OtqFTuenOIKRAe7VtSJYRjM+vkoooIMKBGJlwZZxryyFMg9I2275SlAWz8cJ8WQgCAIgvABFAVEmLKA8LQPRInMAnGjhLVAFJRV4Z7Fwvx9Qrgfqk0MrhWU42pBucsC4sDFfPyw7xIAoE8LIdNiaaUJAYYGdptK3w6Yq4GQZsDwN4D8S9YVN+sJDezMEARB1E/KqqyTQ3HCoUdimKTd41MYlXIBUYlFW89a+SUEGnQAgGsF5XbDTmvi0g0hiqOsShBHJRXVDU9A5Fj+h82S2ZoX9ZgGdmYIgiDqD2WVJkz6ai/6t4pEeKC1CZubwmgS6ocJNyXg+3/Yp3RPCwi5BWLTyWxsOmntKBlk0IKbYSishYDILCznl3NLBH+L4opqRLu8Vx+looh9tZdtsp5ATpQEQRBe4rfDV7E3/Qbe3Xga//nlmNX2SJGomHd3F3z18E0APO8DkVNU6VC/8AA9QvxYK8T/rTgosSQ4w40S4fNyi4XlBZvOKHWv35SzIbow+k5VTVchAUEQBOElSivt17SQOxAaNOwl29M+EJfyWCHw5JDWdvslhPuh0iSImf5vb8WJq+wNMqe4wuEMleL/g9gX5LfDVx0ec72hwiIgfKgst6uQgCAIgvASZXbyORh11pdno55NKlVS4T4BkVlQjozcEkkbZ0kY3jFG6S08TcP8cb1IGuI56sMdeH/TafR6YxNufXcbjlzOr3EMSg6kHHN+O25zW72ELBAEQRBEbckSzf3L4RwUxQQbWbc1eZ4GV2EYBnd/sgsD30nFqcxCft+FljDOpIgAnHlzJMIDlEMMm4T6WWWlBID3LVMP6TkluOOjnci2c5yANAeGnK92XmhYtTbIAkEQBEHUltwS274GnFgQw4mK4opqt9xUK6rNfNXPlf9cxqnMQvx7FVvjwl+vQYBBC51GjaQIf8X3Nw33w+tjO9X4Ofd/scfudnsWCAAocpNg8glIQBAEQRC1RSlyITbYCL1Gjbfv6WK1LcgiKkxmRhLu6CpFooRROcUVGPH+Dvxq8TsQWx1C/W1bIG5KCsePj/Wx+zlns4vtbrdngQCkjpX1HprCIAiCIGpLYbm1gHhmaBscnTMMvZLCrbb56zXgylQUldf+qVw8FfKrzGFRbOAQR0l8+lBPfjnIyFpEmoZJE10pUWWy7e9RKsuBISvFgT3nc2vcv2R/ldV47Jv9+Hn/ZafeVydwYZxkgSAIgiBcRckCEeyng0FrXYETYKMyAi2JldwhIOT5HsRcEZXxvq0LWx2ya0IohnWIwXPD2kiERExQzRko71m822a9DPkUhrzbrFVHkV1k349CzOLUc1h3PBPP/njY4ffUGfwURpB3x+EGSEAQBEF4Cc5ZcdbIdnxb56Yhdt/DPfUXKVgvnMWeCHk6pQ2//ODNifjwvu5YMTUZKpUKM25tjeEdY/ntapHJ4PFBLfH4oJZW+zt8KR+rD12xar90oxTnr5dYta+fOUCyfiHH8RwTxyyhpD5HVTlgslhzaAqDIAiCcBXOAnGzpf5DoEFrVXlTTrAlcVN+LTI/ctiK5vjpsT54fLAgAow6De7oGg9/ve3kxSum3ownBrfE00Pb4P9uVc4fcfBivlXb9jPXFfu2jQ3ChqcFEVFtdjx51lWR9cSn4KwPUAF6skAQBEEQLlBeZeJTUjePDEDqc4Ow7flBNVafjAoyAIBV/gVXsDWF0alJCHQa524PfVpG4Pnh7aDTqOGn1+Dw7GF4QxahoZTbQl75U0ybmCB0bsJaZMqdcBo9lVnEL284nok00bpXKctnXw1BgLr+336pFgZBEIQX4KYPVCq2pgSXEromot0oIJTCI9Uq1uJQW0L8dHjw5kTEBhsx5et9AIASma9DQWkV5v55yuq9d3Zvwi/7W5Jn1RTqySG3Pkz7Zj8A4MK80Y4P3lP8+Tz7WuGjUyxOUv8lEEEQRD2Ei8AINGglPgQ1wQmImpIzOYLS0/+ef6fUer9iBreL5gtuZRdKRY/YJ8JPJFrm3tWZX3ZGQFzOK0XfeVsUt10r8PK0RlU5cD6VXda6Vvbc1yABQRAE4QU4/4dgo2OWB45Yi4/ERRcLV4lRmsKwlXXSVTRqFRbd3wMAsOlkFg5dyue3Xc4TjuHeXk3xQHIzfDmpl8QC4mcREI5MYbzx+0mb205ec+2pv8pkxhu/n8ATyw+g1xsbsT/jhkv7wT9fCMtTlUWOs4z/dDf6zt2M3eecC3N1FyQgCIIgvAAXgRHs4NQFR/eEMADA9jM5tbZCKDlRapywhjiKeHpm7KKd/PIJ0U29ysTgzTs7Y0h7af0NPx07016TBaKi2oR1xzNtbs8rcc3p9Mu/0vHFX+lYe/QacoorMWXZPuQWuzB9lL6NfU2ZA8R0dGks209fxzMrD+H8dTYx19WCMlwtKIde6/5z5ggkIAiCILyAYIFwzhWtY3wwmoaxNSiW1LJOhDtySTiC3L+DYRiYzAyOXCrg2wL0yn4Xjk5hnLpm31FSKWmXI2w8kSVZzyutwv2f20/NbUVFEZCxi11uPsB+XxswDIOJS/Zi1YErWLbrAgAhQ2d4gMGlfdYWEhAEQRBeoMhFC4RarcKUfs0BAIu3ncO8ddZOiI5QXmXCzwfqJlOjXEDMWHEQP++/zDtx9kwMw3SF3BGA41MYZ2pIl13gYthrfql1Gu20LCejOtJ3AJXFQFhzIK6bi+MQxn8lvwzlVSZeVEUEunfayVFIQBAEQbgZs5nB2exiu9aBAhd9IABpsqlPt513yQohf7L2JHKRtPbINbzw8xEAwK3tovHz9L6ICFR+iuacK0sr7VtLnrORdXJoB3ZKxFUBUVNYbY2krQO+v49djmjpcvjmtQJhuup6cSVfiE2nUSHI4J2AShIQBEEQbmbBptNIeW8bPtx81mafvFLO/Oy8gOgYL81W+eexTKdFhLkOS2Tbu8HV5LTpyBRGZbV1kqmP7u+Ov14cjB7NWJ+RwjLXpmts/V/t/b+vF1Wgmqv9seG/wgZ9oEtjAKRRJJdvlPJ+GOEB+tqLHBchAUEQBOFmFm5hhcOCTadt9uEKVIW5EPVg1GnQVWSFeHz5Afx25JpT+xDX4Zg2oAUAoE2M6zc4e9gLU7Xl+8DhyBSGvJpnv1aRGN4xFk3D/PnpE1ctEOVVyhkwl++5qNiellmEm97chP9bcRAwVQG5Z4SNWtd9FbJFeT9ySyp5cZoUEeDyPmsLJZIiCILwAnkWARFuo1R2TSz7V290e20jv/7N7gu4o2u8w++/YYlKuK1LHF4a2Q5dm4aiR2KoS2OpDf41mN+FKQw7AkIkLv5+aQgf6goAwX7s/pUKlzmCLefL/64+hn6tIpEUKb2Bf7b9PADWKoSiaOmbzK6XYM+T+WJsOslOQaXIolbqErJAEARBeIEbpa5bIAAg1F+PAW2i+HVns0feKGGfaBMj/KFSqTC6SxziQmouy+0qvzzeV7HdUQuEvZTXnH9EkEErEQ8AamWBYBjGZr0QAJi16ohV/3PXRc6chdIS6TC7HvVSUKo8/uhg70RgACQgCIIgvIIQgue6B71eVK8iwE6hKyXyLDekUL+68eDvbvFFkGOvQBe7nRUQ+zLysD8jz2r7qcxCzPr5KABBbIjhBIQrYZwllSbIXR0+n9iLXz5+RZqc6tfDVyWJsqou7pW+mXHdAsFFYcTIBEOQk2HA7qROBMSiRYuQlJQEo9GI5ORk7N2712bfpUuXQqVSSf6MxoaR9pMgiMaHUs2K0spqXLJkYazNHDZnRQCAanPNTpFZheVIzynhxwAAgV68AQFAgKEGC4ROGN/zP7GRFvmllXjsm/14a90pTF7yD/ZeYLND+tsREK5YIOSZOif3TUJCuGClKaqo5qeiAOA7mV8Ec2SldIdumMLo0jRU0h7kQhSPu/C4gPjhhx/wzDPPYPbs2Thw4AC6du2K4cOHIzs72+Z7goODce3aNf4vIyPD08MkCILwCM8qhBeeyiwCw7B1Lbjqmq7w2EAhd4JYTCjxzvpTSP7fZgyen4ozWUUoqWBvZko33bqkJguE2KrA3ayX7rqAdccz8UnqOWSKsnH6KeyLC5MtrTShyuR4SfA953Nx5DKb6CpAr8EfT/bHf0a3R7vYYDw3rA3f7+z1YtwoqYTJzCBSci4ZaG+cke7UEOzw58vhyrf3bRkhaW/QFoj33nsPU6dOxcMPP4wOHTpg8eLF8Pf3x5IlS2y+R6VSITY2lv+LifGekwhBEIQz/H1eWpfgoMzsbjIz+HY3+1DUId71GwoADOsYi28fSQYgTEmIMZsZPtxw0dZzfPt3ey+i1OJ46OzUR23g0mSP69WUb1MKwRQjFjic6Mm2UYnUT6FcuDgHhaNWiKv5ZRj/2d+Yaqki6m/QokN8MF/ifMatrdEzkZ2S+f3wVfR4fSMWbDwtsUZEIR/q6jJApQZuex+I7wGkzHbo85XgElq1jg5CpChxlCt5RNyFRwVEZWUl9u/fj5QUobqbWq1GSkoKdu/ebfN9xcXFSExMREJCAsaMGYPjx4/b7FtRUYHCwkLJH0EQhLf4ame6ZD1SZmFYvicDqw6yVSjbx9VOQADCnPiNEqmXfpXJjFEf7sD4z/62CoH8af9lFFlupv41TCG4k3VP9ccLI9ritTGd0KNZKDRqFXo3D7f7HqNWGF+1mRUbtlJwK1kzNGoVX8GUm76pifPXpf2UrDTcTXyZRQx+tPUsfw6aqbLwoHYz2zG4KdDrYWDaViDY8SgZObzPir8OkaKkW960QHj0k3NycmAymawsCDExMTh1Sjn9atu2bbFkyRJ06dIFBQUFmD9/Pvr27Yvjx4+jadOmVv3nzp2LOXPmeGT8BEEQzhIbLPXZkt+8ObM4AAxoHYXawjlhFpRVodpkhtbylHw2uxinMtmUy7vO5UjeU1Rezd+E69IC0TomCK1jggAAy6fcjOKK6hqncMQCx8ywVpUiJx0iezQLw7rjmdifkYebkuwLFgDILpIWKfNTiHCJVMicyQmP7Yanhcbw5k6NVQmGYXgLRFiAXjIVU5fnT47PRWH06dMHEydORLdu3TBw4ECsWrUKUVFR+PTTTxX7v/TSSygoKOD/Ll26VMcjJgiCELhheVJMivAHYO2Ix2UU/Peodugjm892hRA/HbhEhOJpDHHthG//Vk56BNTsxOgp/PQah/w/IgMNePueLvz6jdJK+BdfxP+0XyBRJa2+qRSFAQDNo1hH1cyCcuzPyKtx2uR3WVIupf0qjb1SyccivIXdz3KEkkoTqkzsVFSYv04SWmovSZen8aiAiIyMhEajQVaWNOd6VlYWYmNjHdqHTqdD9+7dcfasckpYg8GA4OBgyR9BEIS3uG55en0gOREAFwooREhcy2e3d24S6pbP02rUfKSBONlQlsi5cMsp207rNTkx+gLjeiUgwmJpyS6swAt5r+J+7RZ8rZsn6Rcfohyxx+WaWLrrAu7+ZBfeWa9sAS+rNCE1LVvy/9KiGtNKPwNOr5f0bRrmb/V+DUwwQFZ8K0K5SJgzcL4Veq0afjoN5tzBlgP/7+j2td53bfCogNDr9ejZsyc2b97Mt5nNZmzevBl9+vRxaB8mkwlHjx5FXFycp4ZJNDTKC4Hl44BDK7w9EqIRwj35Nw1jw/1MZgYVlidehmFw1WKBiA91X3h6qEKoojg6geORftbmdG9HYdRI5jFg7bPoEXAdADu9kGRmLc2JauFGHxmox/RBrRR3IY/O+HyH1E/FbGawfE8GUt7bhslf/SPZNkGzFSOLVwPfjZO0cxYmDi2q8YP+daQZJ0s/vMUge0fnENx3KsxfB5VKhRGd4nB49jBM6V9760Zt8Lj0fOaZZzBp0iT06tULvXv3xvvvv4+SkhI8/PDDAICJEyeiSZMmmDt3LgDgtddew80334xWrVohPz8f77zzDjIyMjBlyhRPD5VoKOxaCJxZz/51u8/boyEaKIXlVQgyaK0KGXG+BdEiX4iSimqsPXINl/JK+doK8oyJtYG7QYrTPSulbpaXfVargAAvVXJ0mMW3AADuD7yGjXhAMQKjX6tIfP2v3jbN+fJsl3Kfhj+OXcN/fjmm+N4WKuUaIy2ihLohahUwQHUEvdTWtU/MUR1r/aTOWZbCRGnP5SXSvYHHvznjx4/H9evX8corryAzMxPdunXDunXreMfKixcvQi0qb5qXl4epU6ciMzMTYWFh6NmzJ3bt2oUOHTp4eqhEQ6HgsrdHQDRwdp3Nwf1f7MHjg1rihRHtJNu4+ekQPx38dBqUVZmQW1IpyQcRFWSAQeu+J38ufFFcVEqpdsTgttF4e10av94mJogPTfRJygWH06YM6+/AnNlo1S0iUG/XF0Duw2CQhXseFTm2ytFBOeIjPECP54a1wcp9l9GnRQS0BzdZ9fmmOgW33CiViA1X4AREqL/3RYOYOvnmzJgxAxkZGaioqMCePXuQnJzMb0tNTcXSpUv59QULFvB9MzMzsXbtWnTv3r0uhkk0FMT55s2OJ44hCEeZ89sJAMDHqWxuhct5pcguLJfUTggyavmn+2NXpDeo+FD31pzg/BjKqoTvPpdpUkxUkAF/PtWfX69tHgqPc114otdrVNCjCuPTnrbqdmf3JnZ3I49UqBI5UZZXmfCppQCWnKQIf+htCAiAzQex/YXB6N8mEiGQhn6aGBVmV0/GyWtFdsfmCMIURt2kHXcUH5aeBCHDbAJyzsAqOb0ck8iJqYCicgj3Iw6jKyitwvAF2zHigx0oLKuGyZJSOtCgRYilEqS4PgJg29nPVbgn7FI7FohHB7ZAZKBB8hQb78HiWW5B9PsNZIrRVmX9e348fB8GNbN/Y5X7eYijJdIybd/gVz1+C0a2V67hISYiwIBgVamkLY1pBjPUSM8ptvEuxxEsECQgCMI1/nwB+KgXcOp3u93Kim4IKx90ASpq/wRAEGLEN6D9F2+gpNKEGyWV6PraBr7dX6/hL/gnrkoT3LnbAsHN6StNYbw+piMOvzIML41kPfbFxbN83v+hUrj5+lUXIlBVZtXlhdL3gNWP292NvGQ4FxIJCFVRlQgP0CNYJ3pg2b9UsV9UkAHBkAqFw51mAQByim3v31G4709YY5zCIIhawzDAP1+wy38tsNktr6QSZzJkTyl/L/bgwIjGiDiPgFwccKhUKj464kKu9Ok0zs0WCO4JWywgLt5gPzPUX48Q0Y3HKJr/12m8l0PAISqEm7K+uhD+sI4sAQCkrQVMtpNLKUWacFakzAIb++QQWzR/e0qxS2KEP0JU7BTGd+FPAC+koyyejTS8Xmy/RklNZBWWY8MJNhVC29igWu3L3ZCAIOoHZaJ6AmG2M7udzylGMKQXa1zY4aFBEY0V8RSGvSdM7sadI7uJcCGe7oKfwrBkvTxxtRBns9mbr/zmKY4aSQi3zmXgU1QKfgXaqmLcobFdAgHFWTY3KQmI7/9hHzTEAsKoUEvDEQumTqPmfSCiomMB/3A+hXmOjbodjpL8PyENwqjOvpXOgAQEUT8oFU1LVNqeUzSZwT8JfFB9F9tYeMWTIyMaIWITuLxA0709m/IFrsTTBWLkJZlrC3eDLLaEkC7fkyHaZj1Nsej+Hnh0YAsMbe/jhQorpTfvMZpdAAAzo2A5Kcq0brOg9D94efUx/Hr4Ku/0+tjAljj52gh0bhIi7VgsS8IlL8m980Pg81vROzgfADC4BztVxNXKkIvH2uBrETO+NRqi/lGcDWx8Bcg9V3Pf2lAmEhAl1212KyytRLDlSWCd6Sa2MfecxBRaX6gp3S7hPcQ+EL8clArUZ4e1Rb/WkQCEQldyPOUD8c3fGag2mSV1GpTSK4/uEoeXRrb3ahpkh7Dxuz0QkmLdWKScrwGwnSxr+d8ZfK0So04NlUqFN+/sBKNOjVnDWgFXDgB5F6Rvkl9/Nr4MXNkPfQn7PdAGRQMQxKOtwl8NARIQRO1Y+wyw8wNg6WjPfo7YAlEoulBUlgLL7wX++ZLtVpwPjYp9OjzPxKFAGwmAAbJPenZ8bub3I1fRcfY6rDlE1hNfg2EYyRSGHHGUw4Tezay2u6MCp5zOIovG1fxyiWjoIn+irk/YsDaqDAp5FexYIAxa6a3uvt4JAIDTWUUoswgIToR1aRqKo68Ox2OFHwKfDwZMMguCnQcYAEAAKyD8FSJjGhokIIjacWEn+2pH/buFlROF5aKrbIKZylLgwNfAmQ2skAFQUZTLvjJaVECHGxr2SRClOfI9+hyV1Wb8evgqvthxHjO+O4gqE4Onvj/k7WERMgrLq+1GEhtFWQ5D/HSYNkBIN3x/cjN8P+1mt49pYJsoNLFYNS7eKOVrJzx8SxLCAnwr9M8pbFggWsRHA4NeAjQGIN6SJ0j8kCFDni109u0doVKxxceu5rORHeLzptOogUPfKu9MPqUhJ4C95nBVREsqqyW1UByFYRhU2xGqvoCPx/AQPo8hSDq94AkYxvopYMV9wKW9QNuRQtuyO9CqnC3SlosQACqUwuLtXilN8uKLjPt0t1W+AML3yLcT9qfE88Pb4oHkZkiMCPDQiFjaxQbhSn4ZLuSW4IZFQNQ2A6LXqVR2YAwLCwcGzQL6PQ2kzgOuHpQ6WteAUadBsFGHgrIqXLUUN1N0oORoPRy4uBuoKAQ2vwa0GsK2yxPVBUQBGtYCxSWvYhigvMpss1KoEpXVZjz4xR6cvKYc4eMrkAWCqB0GUViRp/ItiC8MaovmzdgJmKuAk78K29K3ofu1HwAAxTq2THIJwwkI3/aBqKg22RQPrjy9EJ5DXDLbEXQatcfFAwC0sYT4Hbmcz2cuDPexxENOY0sUBFmqOWsNgJ8l0ZOTDzJcTgUuCsOos3ODbzcK0FkiVq4dEtrl15U7PuIXxfU2ShSygtojNS0bey/cQFGFb/tPkIAgaoc4RvpGuu1+tSHf4lEeEA10vMuhtwSEs+FORYzFmczHnSjtZcP7fMd5t3pyE7Wj1HJRDzZqodeoEWjQ4sRrw/HowBZ89IU36J0UDgD450IenxwpLMC3Eg85TVk++3rPEml7aKKwzAsIxy0QAPipHc5fRCIgqmQJq3T+wOB/W+9ELCDuWQK0HcGvqtUqwQ+iwjE/iL/O5OCtdadwJtu3r1ccNIVB1I7SXGH5+Cogrov7P4Nzjgppwv45gCE0DrgIZJZp2W+5j09hcCZnJf73xynsPpeLrx7uXYcjImzBleZOjAjAe+O6IthPB3+9ls/06C26NGWdJS/klvBPv75WO8FpOFEQ1w0ISRBSW4e5LiC48Er5/0YiIOT+FCo10OEO4Lcn2XVTFTtVwT2YGEOBTndbfZa/XovSSpPDFog5vx2vN+IBIAsEURsYBigXzdGl/emZz+GeQvzCYApyTEAgPAkAUIL6MYVRVoOn9ta0Gjy/iTqDC/szaNVoHROEmGD3ZpV0lYhAA6KCDGAYwfM/3NccKE3VQP5Fx/pWVwq/W78waf6FkARh2REBkXcB/2rF7msWl9JblhbaKI7UED8YBUQDrYcBeoXpWu7VoJwh0tlIDHvi4fOJvRzaR11CAoJwnUt7WD8Ejuunai505SwVRcD2d9hlYyhK297p0NvMrYYDAEphmcIoz3fvuNwMF0omxlbsOuFdZv96HEANc+ZeIilCmlnS18o/4+d/Ae93Bo79DCzux/7Zip7gBYGKfcIXR1JoRMflHy7rL8NUDXzQFa9cnoZfHu6Au3uwDyFyC4TEyZETEFHtgaePA8ZgQKMV/CAqLA9OXJK6gCjFj+YTfLnBl2FoB99L+kUCgnCdJcOt2zbPce9nbHwFuGFJUuUXinKN/Vzwf5h64/0m78GvaWfphgNfs080Pgr3hDKkXTRGdIzFSyPbYfmUZDSPZJ3vdBoVzGZypvQ21wrKkG1JTWz2QedWcbXGQIMWBq2PiZwTa9jXn/4FZB5l/5bfq9yXEwR+oYBazU4jKCG2QMijIgA2QsNC9/BKPqRTXphKOoVhERABkYBWJDQMlhwenOX1ehr7GtVWcWhRlnTW1x1IZ202M1D5eF4vOSQgCMcwVQObXgXObbXfz06hqxq5ngYsu0PILVF8XVr9zhiK8ioT0sxNAQAXdC2FbcnT8Xmn5Xi86ikUxvZBkFGHuBAjMswi1V5sO9FMXfNx6lnc8dFfyC5iPcA5s3iQUYvFD/XEowNbonuzMPz5VH8AbOrkqV/vw950D4fMNjJ+PXwVr6w5xpfgromcIkGEVpt8T0CIoy56JdVchtonuLJPuZ2LquAEwuh32df+z0r7GUPZV8asHPZ5ea+wLHL6lpfGDvETCQrOKsJZNzi4qQpu6uLKfvbVhoCItUxvZRZYVxGVw+aLkLY9OaQ1tGqVT05fACQgCEc5sJQVB9+M9dxnrJwIpG8Dlo5iq97Nb8VeFDiKs1BeZcL0qplYUT0Y/yr9P5TdvwZodxvQ/xkcq24CQIXYEFb1Tx/UEn+YRV7xcs9qL1FYXoW316XhyOUCfLGDjVzhfCDkseIGrRpaS7rhzaeyMe5TO8WECKd5csVBfL07A38cdSwRWmahUHipSulp18uEiqIuuifUEwEBANUKT+i8BcJyE28zHHg2Dbj1ZWk/nVGYWlCaxji3RfFz5P4hUaL033zlX/8I6b54AVHIXk/ObmLXWw+z/lwAcZbkXldrqvgJ5WmOp1Na4+ArQ31y+gIgAUE4StYJYdlO2dxacf2UsCy2PHA0uxnlVWacZ+LxUvVUnDdFY5epPTBhORAYjWuWhDCcU1ugQYsK6JHLZaP0kUgMcfnnPRaLQimfTlcaGKVSqRBgoGApT3AhR/g+2IuCESMWEL5ugahXIZx7Flu3ZbCFs3gLBMDmf1Cy83N95P4UN84LN3lA8hARZJT+rtRqFZvy/s9ZQI5lasKSlprHaJnCqCgC8i+xPmCGYCCmo+JhcRaI7MKaBYRSzQyVSoUgo++eRxIQhGPcOC8sf9xH6k+gCwBaDBbWv5sgRE64i6GvAd0eQHm11NnwmZWH8XHqWTAMg7Qs1qzY0pJ9L9By4y3jIjGqZGW+vYR4PvTo5XwUlVeJLBDWP8nABi4gisqr8Om2c7h0o+7Oz+W8Ugyan8qvz/71OA5erDkM8GKuIDrs1cPwFuKnaolJ3hew5zOSdVy6bjYDuy1JmRxxgLYViSGvgVMt3MgTwgSH07ssjpX49h5gzydC/5umSN8vtkBwDpTB8TaHxYkURwpqXcn3DQupM5CAIGqmOBs4L/J9yD0DZB0DYHkSeHSbNMnK6T/Z9LKO7NdOARwJvf6FokqzVbhjQRk7HXDpRhkKyqqgVavQOoYVEJxyz6uy3IB9ZApDLCDMDHDpRpkgIBQ8+xu6gPhg0xnM/fMU7vx4Z5195ppDV63aZnx3UKGnFHGYXbUPOrXGhQhVPn1OQCj9/pJYHx8Uys6H2F+p0IHpJVsCQh4yKhIQSZEB+PpfvbHo/h54fUwny2ddFnVW8XUteAyWwmR7PhWmc4Nth5Zzv11H8kBsPJFVYx9fgwQEUTPHfrZuy9gJwHIBDW4izEFy5NWQldJsAhYlA++2Zc2BSk8nnMczgA1nitD51Q34aOtZxd2dsOSMbxrmx3ueR1tKKXP1MAoKC+yPqY64LssqmVlYhouWp28/vbVYCDD4mCe9m/nnAmt2zimuRGG5h6bHZCh5xV/JL8PcP0/adag8KxYQPmiBiA8VclL4nIDIPSNdV6mBgS+yy1y2WQ5xVlvOedIefqHsq5WAuCRdr5JOJQxoE4XRXeKUpwl1ftbTJZwFIue00BYUZ3NY3H5LZJkozWYG8/48hV8PC8IpywE/CV+DBARRM3kZ1m2cv4JKzf7Q9DIBUZO/QVke72VtzvhbeXohsg2/+MSKQwBgMwrhfA57YY8WJfXhqhOWM6xZt7zENwrT5MoExGPfHsDu82zYWPdmoVb95Re3hlYbo0mY8NT8w95Ldnq6D7mI4/h023l8seO84rbyKpPEzOyLZyE+VPhf6jQ+dHlPWwd8OkDaFhgLhLARVSiRVcvlpjRa3ipJD20TztFSPnUqL70tThClRKDIWVHpmqSUMEqptLgFTvzLHSS3n7mOxdvO4ckVgtWroKxuxLM78aFvWP2h2mTGz/svOxTb2yBQ8mzOtVxkuQQvOlmxoBoKa5lLhB/yib2bkHtdar4rZoxgRMliqmpwWDuXzQqW6CDBk5qL6y4HKyBCDn9mdx91BXehsARXoLJaeJLt0czac14+hfHt3wqCrh4jdhyti3ngs9lFWHvEtll87p+ncOyKtbUqPaeEN5QFGbR4996unhqiyxh1GtzWJQ7dm4WiXaz9nCl1yorx1m0hTYWph6pSab2atLXsq9i3yh7cfnYtlF575CJg3YuSvBBW1PTgYwy2bpNbX0XwUxgyAZElcqrkrgdy6xtX28SXIQHhAnP/PIVnfzyMp76vec60QaAkILItTwic+U5ugZA/UWx4GVj1KJ/opTRfeDIwXtuDyR9v4NfNjAqPVz0Fs8r23P8trSIQIXIYO3fdYoEIkqYV1mvVuFXNnidj7kkssjEFUpdwDlVdmoY61H9wO6kn+Mtr2P89wzB4/sfDePGnI/XaKlEkunDasgy4k692Xqixz7pj1r456ZaojW4JoTg8exh6+egF/qP7e+CXx2+B1pcsEEqENJFMU+K7cezroe+A86nscrM+ju2LExAVBcCvTwrtSoLg6kHWr+LrMcDZzUJ7dUXNKe+VLBDya58IznpYWmmSJIK7USJ85znnYU5IfPJADzw6oAUW3t/d/lh8AB//hvkmX/7Fzs/tOleDOayhwAmI3tOAGIuzUbnlCS3YIiDkFojCy2wiKIAVDbs+BI58zyd1KSsQBERoSToCGXZ64bw5Fl0qPsd2c1ekXbc9J5gUEYDU5wfxFgdeQAQbJP2CjVpUiWrGvbP+FLzJ0p3p/PemZ6JjcfrjeiVg16xb+XWDJWf/t39n4Mf9l/HDvkvIdTAM0RcRe6jn1IFVL9iGb8CvM27hl8MUakict3zHWkQFsCF/RM1UlbMOh0oExrAZJjkydrLOlKunC20OFs+ThHoeXyUsKwmIsjxgyxusSPlWVN3XVkptMUoCwgELBCB1pDwgivjhrG6cgOjUJAQvjWrvMzVW7EECwkmUTJv1leKKanySeg4ZuQ74KwBA+zuAlFel27gfrlptnUyFq5xXLTJLW8JBK4oF8RWJfKzQvwkAuI5QFIP9QR4sVDAXWgj20yHIqEObGPYHzd2ExFMYABBs1GF61Ux+3R+1v0Gl55S4XF771d+EfBr9WkXa6SlF7Aehsdy8OCELOJ7HwBcprGMLhK17f/u4YEy4iS3SVFpRjbJKkyS09Px19nfChQkTFsxm2yGaR38E/nxBeZtB4ffNWR44Ah1MoORnQ4wr+TGU3oDEg4Wb8riwQ9ZR4YvCRWGIsSMgDFo1/3vlHCkLy6uw6aQwZVtYVoWM3BKUV7HWWVsC1xchAeEkhy7l88tadf2uT/DehtN4a90pjF1UQ/gcnxEuzDqsKaqdsNxyiOx9FkVfKfoRF7BhUtWFyiFLf5iEzJHvVI9DaVIKDvX9yKpfsCVEU+4fIJ/CePueLthm7oJqhv2qB6F2uQau5Jdh+PvbMeL97bX2gYkP9cM//0lxqK9SeGe26PMbjICoAwsE96T34M3N0D5OuInpNGrBa77ShPGf7Ub/t7ciLZO9wVy1pCNuKnL6bPTkXQDmNgXWPqO8vTRHuR0AtAbrNrH1AQDUDkYgyQUEZyHlpiTE2SuvHZZ+dsEVNmnVqqnselA80O9p4DG5oIDTFgiVSoUAWUGto5cLJHqrqLwaP+5jr4utowMRbLQ9detrkIBwEu5iArBx4Ok1Pb37MLvOsT/uvFI73r8MIxUQ3BQGRz/RhUMrM/ue3QJc2gtUif5HnG9EwWUosdvcgV/ORxAujvgK5yOtHam4BC3+shBH+RRGr6Rw9EwMRzHYi36gqqxWom/zySxUVpuRU1wpeYqQs+lEFsYu2slPrSgR7Kfli+3UhE4jPA1FBxnw6q/HJSWC67OAyBfNBxeVV/N1QTz2eZbve/PIQDw+iK2ncm9PNhqAu9iXVFTjyGX2JjT8/e0Y8PZWnLzG/vYjAhw7Z42CvZ+zv+99S9iU0T88CPw4WYjcUkpRzcHdxCd8p7x98H8dH4dcQFz8m33lHl7aDAdGzWeXL+wArovCMIuuSUPVwxJZS2usrCAfoOxEaccHApA6UhaUVeGBL/ZItheVV+O0JQnehN7N+GJf9QESEE5yKU/6BLvzrB2F7ePI6y4oUlkMMJYLul8YW0bXX2SFEM9hamVzdn8vAr4cKvyYAT6rXGGmcp6I80wcRnaK5dfLKk2KOeI5M5+85LV8CgNgy/YWMeyPPBilKLLsL6+kEs+sPITdTviynBIJyDNZtsXBlK/34dClfDz342G+Te7o6Eycvkqlwv/uZC9oF3JLsXTXBcn2+uoDUVlt5s8Hh6etEJwFItRPh9u6xGHTMwPxv7vY/22AjcQ/F2+U8u+T11Bo1Ih9DL65Ezj5G3D8F+CDLmwSp2o7uQ00lt9qu9FA097W2zvdZd1mC7mA4KZPufHpA6R9rh6Q9uVqXwBCSKgSihYI+xapAJGA+GCTkAsjLoS9XhaVV+Gs5UGjTUz9mh4jAeEk3Nw3FyJl7ybi64hvvuJQQgDYd+EGpn69D3tOWEppawzCDyUhGYq0vx0ISbBuFxezKcvD3+dzoS22zgQIANXQonfzcLSOZn9I645nKqaBDeEFhGDu02vVijdlo04tsUDklVSivMqE9zedxqoDV3Df539bvccW+aXCjfpMtjRU9VRmIR77Zr/ESiVODlMh+h+/eWcnfuyRgY490XZuojD/qvA59Yk8y/9To1Yh3nJBddW/xBFMZob3Y0oI94dKpUKr6EA+Z4K/5WJfYMcqFxFIAoKnwE7ejsPf27dAiG/oSjdhW34NNe0LYLPcms2CD4QugPXh4hD7RqS+JX2vvAKnGEUBEWDdJoITEMUV1bxFUq9VY1wv9lpZVF7NOw9zoqK+UH8mW3wE7umoV1IYTmUW4UI9nsJQi0xl+aWVkiRM9yxmqz5ePXkBaw1gf6Bc/9HvspaJ3lOlOzQEAU8dRt7upQjbKJraED+llOVjw/EsTFPZTuoUbNTxKYM/3XYe0we1tOrD3WzEvgHRQQZF85+fToNCi2NmIMr4GgiuxMnni24scvF418e7UFppwjVRjLdJZHUQm+a5iwcAfDmpF178+QheGtXe7mfbu3Fl1GEdCXfCTb2E+esQFWTA1YJy5BR7zppyNrsYeaVVCDJo0UMhaVegZUrMXvXEMH8SEDwVdh6gbpyXCoPEfsDYRcDx1axVsuNYYZuSP4S9G7kcvT9rGeV8LoqzLNlwGUClYacetHp2WiLzqPS9hbLpVHvCRa8UxmlfQIjTWXNWrI/u684XZssvq+StcPXJgRIgC4RTmM0Mf3Hrlch+uS/W0ws3IL0Zih3ZxBnRQlTsBcLsFyo8lQXHAZN+ZS0OctQaTD3UGkfNSUKbWO2X5eFSXilCYfvCEymbhihWsEBw2QvFaZ6Vpi8A1tJSzLD9g1TCWMTTEXILjC3E/iKZheWS/xvnk3AlT4g6EWc7LrMICJ1GJckS2DUhFOtmDsDANlF2Pzsm2Ai9LLb/wZubAUDNkTQ+Sp5FQIT66/mLZ5EH01lzCXyahPkp5kkINLBjuJKn/LsOMmqh19Jlk8de3oSc04IFQh8IjP8GCEsC+s0E7v+enQ7lENfJMIYAz8nSXjvC08eAlDnscnE2cOBrdjmpnyBktA44wNpxioRaDUS0lrbZqYUBCAKiuMLEXy+C/XS8I/jV/HLeqTLYhytvKkG/BCcoKKvi8+S3tTy95onmnssqTdh34YbdXPq+wpX8MompvVB0k+YcegDwN/p/soBur2/A3+dr9hfYd6kIp5mmQoP44lCWh9LSYhhVyjeJ6YNaYkBraaSH3EEw2Kjlzf/i2hHyCAy+PdiIIssURhCUMx3uz8jDLwcv2zVdA0BBqXQsF3NLrcbYKlp4IjEzDCqqTVi09SwOZOQDEDJkOotGrZJEALx8WwdMuIkVEFfz6+cUBndBDRVdUAs9mNKXsyDacl4N87eMwUb1xAjyf5CiJCCmWKYsi7IEH4hb/2vfolAsckh+/hwQGG27ry10fkJU2LmtwM732eUmPYU+SpYOq/3UMI3wr/VA94eE9RosJWIfCO67HeKn431puARlBq3a5WuDt6ApDCfgnPkMWjVvxiypNIFhGKhUKjz9wyGsO85msOuZGIaVj/bhY4B9ibJKE4a+tw2VosdjsZ+B2IktVMV+uQuYADAM8Na6U/jlcSHhji3MjKBN0y9fRXNupTQHeg17sWBUGqjCkoAbFj+Lh37Biy3bQY7cyhMqMiH7i6cwgpUvDvGhRkULhBjOD2JYhxh8NrGXzePiLBB+Og3KqkzILioHwwRjsygiQyWKHzczDJbtuoB31qfxbUohmY4SF2rEecsFJyrIwF+E8ksr+e9hfaKwTDDdBvs5XvrYGXKLK1BQVoUWUYF8nokoG34nSgmkxJADpQz5FEZUeyDIkruhOEsoXlXTjVtclVdTi6fwKEv9HHHuGXG57RocHgHUnP0yIAIIEFkLa/jNcdNixeXVku87Z3UorqfTFwBZIJyCM1EHGLS86dxkZvgEIJx4ANgn2o0nHCxVXcdcyS+ThAAC0qc+TkC0UF1FHzWbNjmfYZ0aqxysQGgWfbWamqRzjG8XsRX4qg2hQOd72Eb/SLZwjgJHZcm7xFUQxY6gtqYwYoKMfHKqQBsWCI4NdkrqVpvM/DREYgS7v38t3YeJS/ZK6lNcKxBPYTA4eDFfsp/aPGWITZyRgXqEWp6Yq82MYrSKr8NZIIKMWr78ursrct776W7c+u42XMgpqdECEV6Df0M4hXBKkVsgxn8DBFisB+YqoMjiLC2P0JKTaHkoiWhVu/GEJlp/VliSsCwWMv5SSyf6P8uGlDa7uebP6T0VCG0GDP5PjV05UfrR1rP8Q1uwUYtImU+Tz1VPdQASEE5Qagnt8tdrECAynU/44A+cf6s/HtBskvR/7NsDPjmdoZQz4HJeGX7afxmF5VWY++dJ6FCNP/Uv4XYN+2SeD1ZAHLtSiOxC2+ZyLseCSfQUrlNJxUoU2LwSjDEE6P8cMPJt4JENkj4xMmuCeO5fHM3gJxEQyhepDvHBqNCy0wq9YqU3b3F+hZooFTlBNg0T5kl3nMnB4cuCyLmQK1g5KqrMfOppfsy1EBBBRi0e0GzCi9oViAzQw0+nQUvtdezQP4Xq3b5RLMwZuKmCYKOOT6DDPaW5g9ziCj6D5KD5qXz2TlsCoqanwHpm4HE/ZzcD300ALu4BTNXWYZqhiayzIhcKyZXTrskCcfsHwIAXgIlrajc+tQZoJUvOJl4X+0DIpx6GvMKGlDpCcDww8ygw0EaWTRFK37VAg9bK2kUCooHDWyD0WqjVQoaxlIKf0KLsCN7ULbF6z6FLCoWo3M2JNcDm12ynk5WhFCb31rpTeO7Hw5i95jjKq8zorDoPg8hPIZ8R5vXFzodyii0ii3Hgq6VRq9mLTfKjQIQ00mL5FGmoaMcmQprhp4cKZb5bRApx07bMz6H+ekwb2g0AYDBJpzA6NwlBl6ZCeKRcuIgprRCcIGNDHHsSrTSZsfqQNGTV6Ej+DRu0rEzDm7olmK79DbGFR6BSqTBH9zUS1NcRtu3f7EW9HsFZvoL9BAtEUYXzFgiGYTBl2T68+NMRSbvc+sNhS0DUNOUo989pVGSfYmtHnP4TWPO41Prw9Ang+fNCMjmj5TfFRUXUZIEIigFu/Y9Q3rs2JIqmWG9+XKr6xEImsg2gtjwIJvWv/efaQB6m/cGEblCpVFbl1utTBkoOEhBOwJVk5bIfcs4xwbDtAb/xRLbnB7ZyIrDjXeAUWwJ359kcnM4qQkFZFSZ8thvf7L4g6Z5rJ87+l4NXAADt1Rcl7e1bNOO/4Pbq1nPz15cZaUSBiVEB476RtKnibJdDbhUtDZfq3CQEc8Z0xG8z+uH+3s349mYR/piZ0hq9k8LRt2WEzf35BbKhWQaT1OTqp9cgPkR4KrH3FFDCW6C0iAq0viDKk1rZHIvO9Z9dVImQQS8wez+QdwH9mP1Ch+veLRbmCJdulOLtdadwNb+Mn+4JMur46ZhchTDOM1lFduvQnLtegk0ns/DDvkuScNnNp5SnpBzNvcFxS6sI7HhhMB68OdGp9zUoLIXwAAA30tksswCg1rFFrwJEvz95vgRNHU79iB0w5SGZYiET0wkY9zXQ+V721UOIxeo9PZtiTDchauPFEYLPF1kgGjhiCwQgCFsThBuHFtInwM120h27BbPIJyH7JM5nF+GBL/Zg2ILtWL4nA3+fv4GX1xyXzMtfF12gbane/onSG2T/7p1wcwv2ApFvV0Cw25aYRkjaN5l7Ah3uwNXpZ3Bf5X/wh/lmqEe+pbQLHs4jHmBDGA1aDTo3DbGqhDgzpQ1WPtZHUnDKCmMo+2KSWk/8dBr0bSVc+LiCN0pwAjJAr8HIzrFW2x++JUmyX1vUZgrDWClYtNQ3zgF/vijtkJMGX+ehL/fg49Rz6DtvC9YfZ38fwUYdPwUlduKtMpnx+5GrGLpgO25b+BcOXlS26ImnCrnw5L/O5OD7f5QTHdlLIb74wR78ctuYIPz+f/3w5aSb+MRTjZZskTgNiAS+u5ddNitcD+SFshyJfnAX4gJccgGhEt3ykm5hpyzu/sK5nBNO0iRUeEAJkl1vxVFVJCAaMlVl0OecAMDw8+5ZheyFTiWq7BYBaYKkM9nFkpu327GkhgYAnNuCpl90xoOajQCAfReEi+3jyw9gyV/puJBTwlsgnhrSGv93qyym2UKQihUZp8wJeKPqAWjaDOWfEO2F2XEWiAro8VLVI3x7qpm1NpSp/bHb3BEvqp8BAu3nPfjxMcEbutbe75YiYJEqqYDQa9V4MDmRz6dgz4GPExf+Bi3axAThu6nJuLmFcOFJjAgQLUtjycVFvxxKIW4DvypRyeEDy4DT66Qdrvu+gBD7iHAE+2n5KBpxkbDHvtmPGd8d5NfPZivnHRA7kHLZLTedzALDAB3irOsX2Mv4N6JTHL/cp2UEOjUJqXfhdW6logg4s0larbK4hgcjgywlc0xH94/LFuIIifju0m1Zx4XlxH51MhxxWW659VacHI6iMBoyy+/FqJ33YJh6H+/78OadbGGpUJVwUYtVWdeU92ihI644FQBc3AV9ZR7e0H0FANiaJkyfHLyYj9d+P4Hpyw/wPhCRgXorRcwRoGb7pJq74gvTaPj7B/AK+Z31aTbrFYjFRQljlCyPW7yb9yNw5IIcK5pacHR6wCaWJwx9RR42PTOQby4qr4ZarcKTQ1ghVVxRbdPx9YDl6Zc7/31bRuLWdoK5NNRPh+eGtYG/XoP3xnXD+pkD8L87O+OryTdh3t1CYZ7a3Iy6hdfg43A9DdUms08679ojyKjjo2iKK6p5h+XNp6RTgPk28nSIRfrnO84ju6icF7N3dIvHnd0Fs/EDyc14fwtbbHl2IJ4b1gbPDW/r/ME0NFbcByy/G8g8UnNfDvEURpcJHn3CtyK8ORt5Ed8diO8h3dbncfZ14IvSOj4e5sGbm0GtAib1SZK0i6fSxGn56wt18h9ctGgRkpKSYDQakZycjL1799rt/+OPP6Jdu3YwGo3o3Lkz/vjjj7oYpn0s6nuiZgOfL/+BZHY+VGx1aKYSLnic6cqj1QVLris2t1NdVPSpPHmtkJ9jjgw02FS9/mC9q8sY9guuUask+Rfm/nFS8X3iGP4yCD+OUhix98IN7M9gBZbRAT+AAJFokDscOY2/ZZqiqgStwoT9ctMSEQEG6LVqMIw0k6QYLpeDOOIiTiRywgL0mHFraxyfMxwd4oPRNjYI9yc3w+B20ZJ+tREQYYxtPwAAYK6fwvD3t2Psop1Wxbt8mWCjFoEGLW/l4pJuyckrVRbjYivFqgNXMP7Tv/nptCCjFvPv7YqF93XH2/d0wRtjOynuQ0yLqEDMuLW1Vbn4RonY8gAIvyV7SJwVaxma6SxaA/B/B4FHNlqLhLaj2ERVg16q0yG9dkcn7P/vUHRNCJW0iy2VTephmXiPC4gffvgBzzzzDGbPno0DBw6ga9euGD58OLKzlZ0Ld+3ahfvuuw+PPPIIDh48iLFjx2Ls2LE4duyYp4fqEAZVFUJFN92mYX6IUQlTBWIBwT3dl1U6ljvBJS7tUWxeqFuo2O6n0/AWiIhAg00LBCcgSiBYEVpECib6E9eUa1mI0xCXigQEt58CS4ieUVvzTVQ836xkhnYKQzDr7AUApYKViDN9a9Qq/vjOXrcdZSInIVy4AHDfC6V5cnF0hyPHbpMy+1E9TP4lnLtegqNXCupVie9gPx1UKhWGd2B9S/6yVLmVfz+VSs9fUkgnn55Twp/bQIMWGrUKt3eNx7heCY3bj8Ed3PqydP32D637FIhyv8jDKusCtVo5IZVKxU5n1vF3QK1WKUaJGbQabHpmIF4a2Q7DOsQovNO38biAeO+99zB16lQ8/PDD6NChAxYvXgx/f38sWWId8ggAH3zwAUaMGIHnn38e7du3x+uvv44ePXrgo48+8vRQHeIm9WnE6oQn1GX/6i2Ztmitvgw/nQb/GdWef9Is85QFoqIY2PEeuyzL8Z6oUk5ixWZPZAVEmL/OZu51P4uz4dQhnbH3P0MAsPPBHE1tqGVxCuBS0RQGZ8ngCsg46gew+dmBWPloHySJxItLqFRsDQ9AyHwJaZ2N5pbPyFCYowcEc+P747vxbeL/gz2zuNhpr9JUi+9DucwC0esRmFQaPF81DQBgqhTi8jMV8nUUV1Rj59kcSTKuuqS0UnkKhhMKcaFCiWPA2nqXr2CBOGDDsZKvO1DP6gv4DKZq4OhP1u2thwrLXe8Dek6y7nPzE2yI5G0LrP0QCAmtogPx6MCW9dLPxqMCorKyEvv370dKiqBA1Wo1UlJSsHv3bsX37N69W9IfAIYPH26zvze46aoQjtgyRI0QUXrk/ro0HJ09FFMHtOC97T02hVFwCagsYkOTXkyXbFLDtvmaiyYJNGrRLi4It7SKQNsYYc4yAgWIuJoKAIiJCOe940P99RjbjU0La8s6zl201SplC8TVfFZ8OfoU3jIqEL2bu2n+NK4b+3r1EH9DT24hiCLOx0OpeBcAfkqgXZzofxWgx+1d4zG0Q4zdHBIG0fHWKlES5zQb3BToMAYY/S6+6LcDG0xs+m2dygQN2PO7bNcFSa0WAJj+7X488MUerNgrDdOtK2wVn+Nu8oGiugFVJjOqTNIvmpIPBJcoSs6xK6yVLLAextf7BH9/DPz8iLStSU9p8ShbhafaDAP+kwn0+pfnxkd4HY/+snJycmAymRATIzXNxMTE4NQp5Xj1zMxMxf6ZmcpP1BUVFaioEBz6Cgttl4l2mcM/SFbj8/YC+5exVSZbDwNgyXMAIMKcC5TlAEExMOrU0MCE5scXAiGj2apw7qTccqxBsVY53rUqMxJVmchgYmHUqdGlSSgiAvX485jwf/TXa2HQarB8Cpu6NWkWm0eij/qEsCNGKn5GdIrF6kNXbc5Fcz4Qjw9qhepcNWBJW1BqsUBcsQgIQy1yIbhMbGfg5K/A9VNYNX0KVh+8gof6CHH9QtU85Rs8JwTF4kelUmHhfc49YblcbdJsFs751C18zQGjXwAqIDxl61GFMmiwct9lHLtSiD+eEpLk7DjDTg38tP8yHpI5dNUFStaduBAj//QVIDoHSsJb6Xt39rqdipAA+TG4yoFl0vX/ZAIavdT8b7KdU6ZWNS2IekG9j8KYO3cuQkJC+L+EhAT3f8hFqfVDo9YAvz0JrJsFLGS9fDOYGFxiLB75Oexd00+vwTD1PnQ6/TGw1MEUqc5QYbmZcDHXsV0km2+x1LH4/f/644dHb0braGloVYCNaYQgP9FcXVCcZBvnSGnLG54TEGEBesy6Xbixcg6VXBieV8x1oZYEVPkXkRDuj/8b0lriGMo9qdoUEJYU2rUdu8v1FCoKAc6yxGX6AxviWikSEAYI50bsq1Imqn/SIkoWZldHZMuid1LaR2Ptk4LAEQsIbrxqFfDbDFZ8y793ReVV2GqJ1Lirh3JZ5dhg2yGbhA3KC4Hcs8K6PpB9SFHLvvvirI9Eo8OjAiIyMhIajQZZWdKY4aysLMTGWifiAYDY2Fin+r/00ksoKCjg/y5dUk4cUxtu6KWfHZhzyKpPFhOOc4yl6pslmY9Rp0GwuPpjqXWIZ63g5sO5m8kjG3CraSG+rWZ9Fpqq2AiNYKMWKpUKIaKbpVGnhtZGZEO4RhSF0HygZBtvYrYxly32fBcnkxE7Y7Kf7wUBEWIRlwXK3xF7FogqUWikIxEkSnz9r94Y1iEGL45wMTSQO99ao6TkcP/WkTBDjUqG/Z+KBQQgWE7ET+/O1ABxFwzD4Isd5yVtkYEGSY4PrnJhSYWJ9x3y02n46Ay5BSIjtxSllSaE+utwWxep2AXY+eWaKmwSCmTJnNblfgxPHgTu/JQN0SQaLR4VEHq9Hj179sTmzZv5NrPZjM2bN6NPH+WSqX369JH0B4CNGzfa7G8wGBAcHCz5czdrMmvOf5+JMFzkLBAWD2SjToMqRnSjvPyPewdmsUAUMH5IeW8bfjmWi/NVEbjAsIInwRIRwjn3iTOd2TPrhmssoqfrfVZhUFxhqMpqZSc8LlFKsFEH6P2BKZuBKZuxeLL0ScWo9YLxi7NAFFyRZvC0IJ5/lyM2p7sqfga0icJnE3sh2tUnYs7/wZJVkyPUX4/PJ/aC2ZIuWFzDBACuFbDOlOIkNuKMm2YzW3L8yOV818blIOuPZ1lNYVTIvkdclteSimqcuMp+v/30Gl4EVFSbJeeHyzvCChHBsmPQqhGg1+ALO6XZCTuU5UvXE6S1aRDeAug6oU5zKRC+h8fP/jPPPIPPP/8cy5Ytw8mTJzF9+nSUlJTg4YcfBgBMnDgRL70kxOQ+9dRTWLduHd59912cOnUKr776Kvbt24cZM2Z4eqg2KWoyAJ9Wj8Yqk20fBl1oE1xjLM5+BWw9CT+dBv4qkcn2u3HAKcdyWhy9XIBnVx5GllLly/IC4MSvfBKpQ1lmnM0uxtM/HAYAZOnYJ+1OKtaxkntiFoef2kv73EQjs2yI0Ftu/PILPwfntMc/VTbtBTTthZtkjpB20057iqA4QKVhU+8e+cFqMzemIgUnSq5kOwCr6poep6oM+HoMsOUNdl3hvAztEAODkXVok1sgMhUEhDjj5vrjmZj963Hc8dFOd49cwt50awtcfKhUTHHnILOwHNOXHwAAqFUqyXTbp9sFKwZ3TCF+OkSILA1z7+qMg68Mq330TmNFHi4c10W5H9Go8fhVfPz48bh+/TpeeeUVZGZmolu3bli3bh3vKHnx4kWoRSq2b9+++O677/Df//4X//73v9G6dWusXr0anTrVnPzFU8SG+OGF6gfQXpWBuzR/KfZJ6d0Vuqsq4OQKoJAVEAF6DbSQCYDv78OFGVcRaNTaLejzwBd/o7C8Ghm5Jfhpel/pxh8nA+e28KtlGulF8mpoD6BQi+bIws/3CXHvIaLaEgEKWc/+d2dnnNzyLUYU/8I2yJ50AUFA2LJA5MoFhIVAgxZ7/z0Evf/HWpfspRL2GBot60FecBFY/RgQ0QpIuEkYox0fCM4CYdCq6z6PwJGVwPlUYd0vVLGbylIoaGibEJwRZbTmRKhYQIiPMS3L8bwXtSG3RBDTnZoEo01MEB4bKK3CyoVzlor8NbKLKiT/8/Mip8mZPxwCYF0emWGE7yrhAuIU+c36Am1G2OxKNF7q5DFwxowZNi0IqampVm333nsv7r33Xg+PynFiLDe7Yti+6RlDYzG8WTxwEkDRNQDAaONhtNWtsOo7aH4qAODCPNuOlVw+hX0ZCjHuIvEAALkRPYFc0XijIgFDOyDrGHoarwJgnx44C0RT1XWEyXPVA7g/uRmwYynAXeeDrP1OuHDEajMDs5mRFLaqqDbxT+8RCvPOYlHhtbzvoQmsgACAi7skAoKrynkmuxh5JZWSG1JFteMpuN1OtczTXUHYAeCz/70wJAkfp+Wit+okntX9iIyrc4DuTSQCoqi8GtlF5Th2pUASYlpeZfLIMRZXVGOLKC31sA6xfPpwMUpFrrgiZYsf7IHHvj2Ay5ZMoVUmM28JO5NVJLFSlHoy+2tjgJvCuGkKMPpdrw6F8F1IojsA58V9g7HjX6HzFy7sllC7tlumKnbtoTqt2O4K+0OG4j+npali7+gaz9a6B4AV4/n2xIgAdFOdxV+Gp/Ba0WzrnZXlC0Vy+swAuj1g1UX8VFdpSUZUUW3C/604iM8tpmWNWqVYWU7stNlNltK1zhCX+pUls2gTE4jECH9UVputMm1yUxiuOlDWCnklQ4UpDLafReBWl+Penk2x0vA6ktWn0G//TPR4fSN+EFWmzCmuwKgPduBfS/dJKsZ+vv28fK9u4Y6Ff0mmhmzlEfHXayXZJ4MMWrw0sj0AoGkYO0VzOY/1o/hpv5DtMLekUmKlMHkpUVaDgZvCkFezJAgRJCAcgHtyLoEfFrX+Ahg1H+j1iDTEsUlPoYBMhX2T8CrDqwBsTwMA7IXTET6/bl3lrn1cMNBUeLJGJZtoR69VY27LowCAliUHrd6HYz+zr6HNgOFvAlprK4JeJAIqLDfVVQeu4LfDVzF/AyuMIgP1ViW3OdbN7I+vJt/EjtEbiBPfmKS+AiqVSjjXsmkMzqRem1LcLsPIvic2pjB4oVFdgbfvEeasI03ZuFFSif0ia1Z+aRVyLDVRxFauL3dKE5K5i/M50mRPJjt1OsRhl88Ma8OL1maWugE5xZUoKK3CS6uO8v2eG8ZGtjw6oAVaRAbgrp5N3Tb2RokNh12CEEMZVhxA/DR92a8t0NtycR75NqvUK4vYpD6c2c9UAfz1vt19asHGuduap9U6GGZ3hrGOfY8INADJjwHrLc6pBVeAKNYi0b51G4B7cKuulIqE4xbfh4532vw8nUYFlYp9gqwwmQDorIpPjepsHU7H0S42GO1ivSQeAOEpHQBKc602c74h4jl4QAgfFOeNqDMqZZkWHbBAiJ/GGSeeEzyRM6FKwRpgr9BXbIgRZyz5QsS+MsFGHeJDjLhaUI5TmVIL0b/6NQcAvDSqPV4a1d4dw26cHFrBOhlftNTYsSVWCQJkgXAI8U1eEn2g0QKBUWxIEyAtYbtJYYpARJwqF6VV0qfcskoTfj18FWeziyXVJ389fBXz16dJEgFx8MmrRAToNWx4FZeHYe0zooMROVzmXZC+kcuPYMdhSqVS8VaIymozKqvN+GjrWUmf1tFBSm/1DVoNEZaLrbObcjU6rASExTlUybfD41TKMi3aeirkckMUXJZE+zhTk1PsJ+EulLJH2is3HiMSMTEyQdPKknJ9v6z+hcaGxYtwgnyLc/Gv/yf4CdEUBmEHEhBOYm/aAWoNm7FNgamVz0jWA1EuicUHgG/+voAnVxzEyA+2S9qfXHEQH209i+RXfkbaB3dIxwNrXwP+6ZP78YvL8VaJ4vAzLBElVeWslcISfooQ++ZfcSgnl5pajDwCw6doO0qwsIgrBloI4AWEVNzdsNwEvZKUyEpA2LBAhCWxr+v/DXx/H99slv3MHx3YwuZHXS+qgNnOzd0VuOyRYaIoIHtTGOLiV0kR0ggjzhq474IgIMb1oukKt1CUZd1GUxiEHUhAOIldAQFIrRAWrrZ/BBvNvZDLCNvaqC4hfONMIPMYOxdfVS7yLmesUv4CwCu6b9A2bxu/Pr3yKas+/xGbb4e9zr6qRSKjUiQgruxnX39/GljQgTVdqtRW6autDtEiIIrKq7HrXI7Vdp8WECoV0PdJdjnfOiOlP59MStkC4ZVjq5AJCFtm5ah2is0M2LBJjoFtoqz6dG7CipJqM4PVh664MkqbLNjI+sYEGrUY0o61mE24qZnN/mUiy5xcsHGZKrmIjj4tIvC/Ozu7dbyNliqFomQ0hUHYgXwgHCTMX4e80ioMams9ZSBBK5tD1vohfvx7uvplFgAAOChJREFUeKvFRYz8eR72Gp8AAHyg/xg4A+DCH2x9+opCVDf/2e6uB6kPSdZPMdKL8I4XBiMhXOQkyGWPY0xs5kW1WmqBKGTDTXH4O6GNMddYBIcL+5uybB9yiq2Fjk8LCAAItRTQKs5kQyRFUQ7+FidJ+fTSjRLuKdoNx3YjHTj6I9B7qmMm4grpfL/Np8Kk/orNRr0OnZuE8tUp28RYi1zxOXtm5WHc1cN9T/UXLNknyyrN+HxiLxRXVtstsX1ruxis2HsJ0Qohnf6y/CV39mhiMyU74SQl1g8DZIEg7EECwkHWzRyAAxl5GNZRuSYHT5sRwJ5PhHV/tlz0re1i8CLCsNvUAX00omqXVSVAPqv8/UsvAQpTEgCgghmRKumNhKtwOe+uzmgVHSgVDwDgZ8n+yJhZr2r/cKlDXv5F6xwDDsBNYSiJBwCIspMgyyfwD2etMuYqoOS6ZMqGs0CUyi0QpZwFwg35K769C7hxnq03MO7rmvsXZ0vXbU1hxHRQbNYw1TCaBeGo5MeRFOGPbVat7uGKJezykwd7QK1W2RUPAFtg66uHb0JHhUgdeQZTn/+u1SfktXr8I4AAa2sVQXCQdHeQmGAjRnaOq9lZa9jrwITvhIt88jQAbIKcAy8PRTH8bL61wk7yG66yphiuwmXv5uHolRRutR1aveCTwcV1iy0QuWeAX5+0czDK6O088em1aknGS59EpRIujCXXJZv8LT4Q8mJhXIZN91ggLLkWTqxxrL8lMRmPPbNy+9ut26rL0UIUtqtSqZAUIRWbL46UTn/Yi5JwhvzSSj4pWqd4G8JHhkqlwuC20Yo1Q7gpDA6lxFOEi4ijkrR+wPRdiqHcBMFBFgh3o9EB7UYD07YBl/YCXcbxm8ID9AgKCQWKld9qrioFbGS7/FY/16qtGH5oFR1o5WgmwT+CdcIrzgYiWkp9IADgyPfS9R6TbO/LgkEhmdLkvkloGuaH0QoVEX2SgEig6KqV2ZbL1pldKFhXvvk7A4cv5QPwwvQMwwhTTRwBdqbRbv+QnaIJjGYtT6f+AC7vxa0Jarx8EujbkrWIrXmiHyYu2YPDlwugVavgr9fCqFPzCbMqqs1uyUjJWalC/HR8hEttkE9hSNLBXz3ElpyOcrHaaWNHLCCS+ilmoiUIMSQgPEV4c/ZPxs3tkoB9Nt5TWQLA2pKghXLp7EHtYjHnjo42kzYBYJNC5WcAX41gc0NwJaHltBoK9JwMtBxse18WlLJMRgUZMKW/be9+n0PJAmE2Y1DOCtyk0uLI5U6oNpmx98INvLxaKG1c51EYlcXWzm06O7ka/MPZJGAcWSeAy3vRxFCO3S/dighLxcoQfx1WPX4Lfj9yFT2asX4Yvzx+C0Z+wEbsFFdUu0VAFJSx312l74xNrhxgQ5AjW1ltkleR5QVdeQHwmaX0/Mu5bIg14RxiAXHbAu+Ng6g30K+srrHjNKeuKlVsj4LyTX/Gra2s/R7kBMYIy3sWs8WkAKDzvawjH0dMB6D9bfb3ZUEpmVKwsZ59lXgBIbJAnFyD2L3/w48GIKn8Ozy+/AA2nJCGtoV7KpHUpX9YYRAriyiQl1V2Fu77Vp6PuBDp9JlGrcKYbkIisvZxwfDXa1BaaWJ9QJQjkp2Cq/oZ7Ofg9yPvAvC5RcTOzmenm0T4i6wY/noN649TXQlcOyx0yjlt0x+EsEF5AXB8Fbt81+dszRiCqAHygahrQqwzR3JUlSqnwI5Vsc5NV5lI3nESQI3OaACE3AAclkqhfCQCRzNZxU97u1TwcfBacSxXCYhkX8UWiFxpQiy5eOjbMgKhnvDvKC8AvkwBFvdjb4aSbfnCcmRbYOxi5/bNCQh5eWYbcFMEch8QVym0JKZy6LsKAJdF5rki60RfEaIpC96qsW4WsEzk+5F5FISTHBA58/or+FMRhAIkIOqaYNvhcSZ5wiAL0ap8AEBAZFNUQXgCc+iprvdU1togRhcA9J4G+Ftuoj0nA62H1bwvC0oWiHqXCZAXECILhNr2/1OvVWP5lGTrUt6maiD3nO3qUDXBMNIEPmc3SrdzU04RrYEZe4Fu98EpOIdLBwVEgMVJUV4LxFU+2nLWsl8HLRA5Z0TLaVabm4QKVhS+6Na+L6WdShXCEQn7mEUO3Do7PlUEIYIERF0THG9zkz+kYZEzBrNzwLF+7MU8IDgMVaJZJ4fM6UGxwN1fCCGdABAQwdbueO4M8GoBcPsHbI4IB1GyQCjVO/BplHwgRAJCJ/M7CfPXWYsHAPjhQWBhD2Db266No7oC+OlfwvrPsgqu3BSGrdDNmnDSAsHVAnlnfRoKSmuX1rqi2sTXtDigVJZeCfHNX55qHayvTQDKMFGzHoGVNoSCvHYIUTM60fRWfHfvjYOoV9SziesGQExHoM8MzNlegNm6bySbAsBmotzz7yGIDjKAYYDOTUMwoPAqsB7QGgJgVgkWCKcS6BiDgTJLnLfWcrFwQjSIUQqd6908wqV9eQ1OQGSfFNpE/9sglOIGhDwEik6ADAOc/pNdPvU7MOjFmj9Xbqm4uAvIEpnczSLhsuVN4Mx6dtnVjIBKvh526NwkBCeuFWJP+g38d80xLLzPtZuJyczg9oV/8et9W0U69kbxzV8htbJGrcJs7dcYp92GoyXbAOYuhX3YCHMibFNlSUnf9X77TroEIYIsEHWNSgUMfxNfmUZabQpVlUCtYpPjqFQqqNUqDO8YCz9Y5sV1ftjfik1fvdFvlHOfy2WlBGpt4m0aJjhuPjqwBf56cbDEtFwv4CxBhZfZ8D8AqBbqegSppA6tWiWxtX2+sGyjBooV8sRd51Ol66YKNnV1/iVg+9uCc6CrFgjOibZYoc6BAmO7Cz46vx2+6tpnAth1Lgens4Qb+cu3OVghU3zzVyh2BgD3aNlIkc7qC8B346w7kAXCeTgBoatnv2PCq5CA8BIp7a1j+aNU+Qj111uHZXLRGTo/DL3vKWxIWY8ejy9x7gOH/09YruUFtmmYcJHp1jRUIijqDTGdWL8CQKgJIrJGPH5zlMSvQ9FY89d7wnJ1uWOfK7+Rc+JFTMFl6/TVfi46tnECoiSH9deogZZRwvy3wUapeUcQG1pGdIxFdJCDT7U1WCAAQA3RdNmZDfb3QTiG6BpDEI5CAsJLfD6xl1VbFPIV/QuEpwN/aDVqDOt3MyKCnPyhB4hMyI7e7GwgToXcqYmLT8beRqUSSnvnXQB2fSQJax3f3ojjc4bz62ol/wdxJEv+RTYUsyZnykLZU326QgLpk79aJ/wKirHu5wj+4ZapGcYq66YS4umpyFqkiRaXvW8S5sR3VXzzd9BqYr0PmsJwGtE1hiAchQSEl1CpVMCdnwFhzfG53yMAgGaqbDQzWpfHRpXlhl/bp4NQS/GtJtbixRlUKhV+/79+WD4lueY8FL4MF+J6/Bdgw3+k266fkiRSUnSgFOftKM1hQzEzdtr/TC6MVs7D64BRlimRi7utk0cFuigg1BqhuqpC+XI5KpUKL4xgMzkG1SK3hziKY/qglo6/USwgbCU9U2LQS+zvSb4PwjF4AUH+D4TjkIDwJl3HA08dwiE/1j+hmfo6vro+wbqoDXej0tZSQDy0Guj1CHCPk9MfCnRqEoJbHHWM81Xa38G+FliX9ZbnElCMUpWfJwA4vc725x3/Bfj5EeVtka2Fctx5F4CjP9nej7NwGVEPOlC4C0A/y3ktKHM9CoNLYT2sQ4xzlgyx9UCcA6MsD8jYBZgUxhQQDQyaBegt0y/y8udEzfBTGPX4gYCoc0hA+AA5+gScNItKc1/aK+3gLgeniJbAbe8BYYk1920MhDQBom1kLDy/FTCbcUdX1tny8UGytMrVlUClQuIve+W5f5wsLN8yU7rNP0KwEN04DxyURuigaW/b+60JLvvoga+Bk7/X2J2LOHFVQFRWm/HGWtafROesH4XcAsFNCf02E/hqJPDbU9bvMQRJXx2YqmmUlN4AVj0KnNsqbc85K2ShJB8IwglIQPgABr0WuUyQqEU2j04OTp4j2EZm0JLrQMl1vDeuK3a8MBhDO8imEMoUrA+A42b3EFlCMZWKHYtKVn8itgvwyCYgWlot0ynEicS2/s92Pwtc0qfSShPMZucSZJVXmTBxyR5+PbPASX8bsfXAVCmI5xOr2ddDyxXeZBljXFdApQby0h2arml0/PUeWzzvm7HS9u9E3w+yQBBOQALCB9Br1DCLT4X8JkQOTp5DntgrsZ8QkllZDK1GLfh57HgP+Lgv+yTHTV/4RwADRfkfHMy3oJhQTKO1FhaRrYGEmxzbpy1apwAvpLM31+zjNd5cxfUmyqttl5hX4qf9l/H3eUFcOSUgKookobTsAPJrjh7hfIT8QoVaIlxkDSFgI6qFLy8PUBprwilIQPgAVuWx5TchzqxLFgj3I3ZObDkEmPy7YAqXe/NvnsPegP/+WKhc6B8BDP43W0YbsD53Zzaxvg9y5DVKOKpkN1B3iUb/cCFsNfec3a5GrSAgSiudExBrj0hLjw9uF+X4m7kbnD6I/b8CwD9fKGaklNDyVmE5phP7unIicG6L45/dGBBHYsm/Zxzcd4QgHIAEhA9gkFsgxImets9nsxUCrmcjJGwjfuLyj2CnEjhnvNMblMMyK4qFKQwuP4ON8uBYfjfr+1CcLd1HVHuh5oA4x4O8IqrejXUJAqOtx6iAWq3irRClFdYCIre4QjHNdX5pJfZbUlavfLQPXh/TES+OcGLqpcgiPoJihOPe8S6wroYMn8PfEJbFaZhPr3f8sxsD4lov8nBiDrkFjCDsQALCBzDo1FhYPVZo4J5izSZgy+tCuzG0LofVOPAXpeDmxESFxTly6xvA5X/YZbGQYMxSCwQgPN2JxZ/YybJclBgq+TE2M9XENeyUyUMiC0X/54CeDyuPr7YoVSC1AS8gqqTTB4XlVRi6YDvu/HgnTCL/iGsFZej22kZUmsxoFxuEm5LC8FCfJAQ5WoUTEEJcA2OlEUdnN9l+T2gzqeNq9weFMOWc045/dmNAbFHjEpWJ8430fJgN+yUIByEB4QMEGrTYz7TF/CqLM1NpLrDtHWC9LDcBWSDcj9gCwd2IxAmMuPlhcXQAYwaKLTdhf8t7uBt9Sa7Qr0IkIMRiYuhr7GvCTcDDa4H4bsK2kCbAkFdE43OngHDMAgEIZb3lUxj7M/Jwo6QS53NKcPSK4Kvz/kahima/VpHKeTNqgosOiO/m+HSd3D9C5wcMs1gkZOXZGz3i7yO3zKVSD2kG3LagzodE1G+omJYPwIXNnWEsEQFXDwFpf1h3JAuE+xFPHyili+YutOJqllVlQtZKbs6dm8KoKmGf6vT+UqvDKcv59I8EtDXkRRDX1TAE2e7nLErTLDawNYVx4qpwTMeuFKBbQijbr0ro55TVQUzWcfa1+QDgygH7fcOas9EWA5+33hYUy7466tDaGKiulGRa5b/XXDG4tiPZ6TuCcAISED5AiKUs9w3GUv2xyMb8pKsFlQjbiJ0ZI1tZb+dMvWIBcehbYbmtpSiaIQjQ6NnQw9JcVkCIn/guWUIbO91d85i0ojLthmDb/ZyFm8IorllA+HEColL6hH/0smB1KCoXtplFUzyBrmaw5P7XfuHKGRHjurHVbJv1AbpOAHLOANEKRbo4q01VqSDmGjtcUTYO7rvJ5Zzh0roThBOQgPABOAvEDdh52vQLo/lJT+AfDkzfxd70k/qzbZFtgZw0dpm70IrFAIchWBAgKhUbMSHOXSB+D+cgGJrg2Lja3wFcTwNaDHLmaOzjhAUiwDKFUSISEKcyC7HuuFAhs6icdaRMzymRRF8EGVwVEJb/lyEI0ChYaaI7AGM/FtZjbCQBUxJzjR15cTbuf11kOZ+hlFyOcB4SED4Alyb5OmPHwnDX53UzmMZITEfp+gMrgQ+6ssvchbZKVtwKAMbJUkNrLU/N1eXs3Pzqx4Rt3IXaUZ+G8d+wjpvuNCs7GIUBANGWolrXRHkcfjkgrePBWSA+2iL1NbAKS3YEhpEKCFOldR+DgyXTVSr2/1x0jRUQjoq2how8t0xFIVtanksXzn03CMIJyInSB9BaakUXwkbIXqe7gdZD63BEjZywJMERz56AkDu1cr4N1RXAsZ+lN2ruCdAZp0h3z0nzURgW3wCGYUNNFeCSZ126wR73B5vO4NPt5yV9OAvE5TyF/42zZOwCzJbQUEOQ9RMzIPUNqQnu/1xKfhAArAVEeaEQWqzW2U/BThA2IAHhA9zaLhp9WkRgurzeAoc7cwEQjsE5L3ICQl5eG7D2TxBbIIquWfcHWCdKb8FFYVSVsDeQTa8C8xIUE0s1swiI9JwSlFeZsGCTEBIZYPGP4CwQ6TnS6pc1VTRXZOkoYVkfqJwSnMsy6QhcivK8DBcG0wCRC7K8C4KACIwhB0rCJUhA+AB6rRorpt1sO+mOM09ehHvg/uf2LBBWAkJkgbB1QfZmqmBDoCBgruwDdr7P5gbg6kyI6JrATqftu5CH01lS/4+2say4KiqvRpXJjOyiCsn2Pi1rGXqqVisLiI53Or6PSEtGxZwz9vs1FricGJFt2NcTq4EbFuFI0xeEi5CA8DVGzGOd3cQhm2SBqHs4ccA9uSkKCJnTK2eBMFUIiabkuDOvgyuEt2Bf93wqtKmtXaFaRQehQ1wwqs0Mlu68INnWOpo97owbJXzZbo1aheNzhmPvv4cgJlghgsJZotoKy9EdgHu+cu4pmXv/iTVsCGNj5txW4KAlcqhZH6F91VT2NTDG+j0E4QAkIHyNm6cDz50BmvYS2khA1D2cOLh2GPhssHJSInmoIW+BKLdOXQ2wc83uzOvgChGWabLT64S2UuXKosM6sjeWVQelzpPDO8XAT6dBVmEFdp1lhVJkoB4BBi2iXREP19OEZc7KMGYR0O1BNkLm8d1Ap7uc22eHMWwkRtFVILcRWyHMZmktlh6TrPsEkYAgXIMEhC+iUkkz8dEURt0jvtFfPQDsX1rze3gfiAplAREQ6f255riu1m02SpNHBymLgVZRQWgRxYraw5fzAQCRgTUkx7LHot7CcrcH2NfQZsDYRdYRMo5iDBGsELbqPjQG/nwBOLCMXe73DNC0p3UfskAQLkICwlfRiawOJCDqHlcsBUoWCJUod4e3py8AabEpDhsWCC4/iZy4UCOiLGGexy2ZKbl1p5F7XLqzZD3nSFl4xX6/hsw/ovBv7vs3ea20D/lAEC5CAsJXEf+oaQqj7nFJQIgsECUWASGubuhNB0oOpUiGsnzFrqH+1gKiW0IodBo1oiwWB676ZlKEi99Rua+IO5M+BcezrzlnXAwNaWBwYbxNe0vbyQJBuAgJCF8lKE5YphTWdU9NAiK8pXUbJyCuHRYKcokFBPdE7E30/lLrFiCt0ihCboGY3DcJ30+7GQAQKbM4tI5x0UpWcFm6Lh9bbeAckXd/BKx5wn37ra9wEThavdSqSQKCcBGPCogbN27ggQceQHBwMEJDQ/HII4+guFj5YsUxaNAgqFQqyd9jjz1m9z0NEvENLNpGyl7Cc2hsFIRKfgzoeBdw/0rrbdwUxuEVQptYCPqCgACsxVFliWK3YFlRrOaRATDq2CmZWJmzZK9EF60r4sqngHstEGI/okPL3bff+orY/0VcOI6mMAgX8Wgq6wceeADXrl3Dxo0bUVVVhYcffhjTpk3Dd999Z/d9U6dOxWuvvcav+/s3wlz2IaKbTWCU98bRmLnlKTZCIDBGcERrehPQ+R7l/lqZ02Gnu6WRGiE+IiDk0SNyAbFrIXDjPEIHvylpDvYTLhecEyUAtI0J4nNDOI3VFIYb/X3kJcGryhwvE95QCG4KFF4GBs6SXkc0okt/AAkIwjU8JiBOnjyJdevW4Z9//kGvXmxI4sKFCzFq1CjMnz8f8fHxNt/r7++P2NhYTw2tftBiMDD8f8pe80TdMNQiYlPnCW32bkDyp+foDtIIAF+xQGhlxyAWEAwDbPgvACA4rDkeuvlWfPM3m81RPKXRKlq40UcEiqqHOotYQAx43jo9eG2QC7oNLwOj57tv//WBaktht45jpe0mUZVVKjZGuIjHpjB2796N0NBQXjwAQEpKCtRqNfbs2WP3vcuXL0dkZCQ6deqEl156CaWlbsi1X99QqYA+TwBJ/bw9EkJs7rUXJSB3djWGAEEiIewrAmLIK+xr62Hsa2Wx4GTIVRIFgJw0DGgjPLU2jxREg3gKo7hCWvLbKTgBkTwduPW/ru9HCbnYE0ckNBY4cSj/3nJ1RwiiFnjMApGZmYnoaKlpTKvVIjw8HJmZmTbeBdx///1ITExEfHw8jhw5ghdffBFpaWlYtWqVYv+KigpUVAipdAsLFYrwEERtED8V24uIkZvfubLSHL4yhdFuFPDUEVbgvJUIMCY2ckRnFKozAoDZjJ6JYQjz12FAmyg0jxSOXSXKZ5FfWoubEScgPBGhIrdAxDhRS6MhYDaxIcWA9fdWqdopQTiJ0wJi1qxZeOutt+z2OXnypMsDmjZtGr/cuXNnxMXFYciQITh37hxatrT2fJ87dy7mzJnj8ucRRI2I04rbncKQXaQNwYBeVO1SvB9vE5bI3mA4KkssAkJUg6KyCOEBeuz/r3Il2M5NQnD0SgFGdY5T3O4QXA4KTwgI+bkqtv3g0iC5fkpYVvpu2kq3ThAO4rSAePbZZzF58mS7fVq0aIHY2FhkZ0uz8VVXV+PGjRtO+TckJycDAM6ePasoIF566SU888wz/HphYSESEhIc3j9B1Ii41LEzUxiGIKDZzUDbUUBcN+9noZSj1rDHU1XKTmMEREgFhKWQmFqtPO6vHr4JW05l4/Yutv2ZaoQXEB5IsiW3QJRcB0xVtiNsGhqf9BWW5f+Lu78AfnoYGPp63Y6JaFA4LSCioqIQFVVzVECfPn2Qn5+P/fv3o2dPNn3qli1bYDabeVHgCIcOHQIAxMUpP+UYDAYYDLVIo0sQNeHqFEZwPHuzum+Fcn9fQB/ICgiu6qhEQNgPuY4MNGBcr1qKdX4KwwMCQknsVZa411HTV5EnzpKL16a9gJlH6248RIPEY06U7du3x4gRIzB16lTs3bsXO3fuxIwZMzBhwgQ+AuPKlSto164d9u7dCwA4d+4cXn/9dezfvx8XLlzAr7/+iokTJ2LAgAHo0qWLp4ZKEPYR34i0dsSqWEAM/i8QoZBsytfgkpRxwkGclbKiyKq7W2EYoDSHXfaIgFCo5aFUVbUhUiHyBRv+P++Ng2jQeDQPxPLlyzFjxgwMGTIEarUad999Nz788EN+e1VVFdLS0vgoC71ej02bNuH9999HSUkJEhIScPfdd+O//3WzdzZBOENQHFtDQqV23I/h5nqS/EwuIMSJnTwtIH6c5FkLhEZB7FU2EgEh9m+4+XHvjYNo0HhUQISHh9tNGpWUlARGZGpLSEjAtm3bPDkkgnAetRqYsoVdtufHIKlfUk8KoHEC4sx61qxdJHI09LSAOLFGWPbzgBNlgGiq1T+StXZUKWfdbHCUWARESDPf870hGgweFRAE0WBQOzDbF94cGPc1m9mvvly0OX+A/UuBs5uBgkvCtsoiwGx27NidxSTLHaGtRTIqWwRGAf9azzrBrphgERBlNb+vIcBPDflAATeiwUICgiDcSYcx3h6Bc4gLtYnFA0dViWuVSWuioo7ytTRji3/xRbps1P1ocOScYV/DEr07DqJBQ9U4CaIxU1OlV09NY3h6ekQOl665MThRms3AptnscnRH746FaNCQgCCIxozcKVRrBF5IF9q/G2c93eAO6soCwcH5Wu39rG4/1xtkHgYYSwKzxL72+xJELSABQRCNGbkFIq4rO2/OTVtkHgXOb3X/54otELd/4P79y7nMhoojfXvNffMvAmtmANmuZ9T1KlnH2dfwlkDz/t4dC9GgIQFBEI0ZeVKlkKbsq9jZ0BMOoeUWC0RkG6DnZPfvX07Huxzvu2oacPAbYMkI620MA+SeY6cJfJXraexrqxTvjoNo8JCAIIjGjNwCwYVTluUJbeKaGe4i/yL7Gpbk/n0rMeodYbmmXBCX/2FfxYXFOPZ/BSzsAWx9w21DczucdccTuTUIQgQJCIJozMgFBGeB6Hin0OaJyAVuSqHpTe7ftxL+EUJlVC7E0RZqG8Fp5QXA70+zyzvedd/Y3A1nPVLKxEkQboQEBEE0ZsTlxmM7AzdNYZdHzBXaPRG5kHuOfY2poygBlYpNJgUAJS4KiE2vStcv/QN8PRbIOlHb0bkX7nzZK/xGEG6ABARBNGZCmwnLj+4ADJYMmoHRQk4LscnfVAV8PgT44aHafS6Xajmg5sJ8boMz6XMVQG2h0gjLnw8BTq1llzOPSft9mcI6mP46w31jdAfV5eyrvdLzBOEGKJEUQTRmjCHAk4fYImFyZ0ku+ZI4/fNng4Gso8AVOJ6l0lQF7P0cSOzD1hQBPFsDwxbGYPa1osB+P60eqLAsX9kHfH8/8GqBsk8EIK074QtwUxjyEt4E4WZIQBBEYye8uXI7l3ypshQ4vQH4fSZQeEXYXl1mv7w5x/p/s/kXIloD/7cPqCoHKi2lwutUQHCFw+zkoKgoti0IbqQrt9eVI6ij8D4QNIVBeBYSEARBKKMTZW/87l7r7VUOCAizWUjelHsGKM5mLRIA62tQUyZMd2LgLBB2BMS1Q0ISJjEHvgHMVcrv8bXCabyAoCkMwrOQDwRBEMpw4sCWz4A4OuP4L8A3d1n3zc+Qrq+cBBRdY5fruugYN4XBWSCunwbyZfU/uPBSOVz0hRKnfgcKLtd+fO6imgQEUTeQgCAIQhnOAnH4O+Xt3JMuwwA/TgbObQZ2L5L2KZMJiou7hERHka3cNlSH4CwQ/3wBfNIPWHQTsKi3EBHCMMDq6exyp7ul71VbHCt7TAQSktkKn2I+Hei5cTsLWSCIOoKmMAiCUEZfwxz63k+B0e9JrQxcBABHaR6s2Pwa+xrZtnbjcxbeApEvOERWlQJnNgIRLYGrB4W+AdHS93LHNewNYdrl/Dbg6zvY5ZpyS9Ql5ANB1BFkgSAIQhldDf4N+5YA2SeAwqtC24109kn+u/HAV6OVHRKLM9nXmA7uG6sj2HLY5G7+JpGPg5KfhEojWDEAIeSVw2TDR6Iuqa4QHFTJAkF4GBIQBEEoU5MFAmCdIouzhfW0tUDWMeD0OiDjL2CLJeWzUkghF9JZV4hzXgBASAL7WnKdfa0UFfgKSwJmHpX29wuT+mzIc1gUyPwp6prKEuCNaMBcDQTGAEHx3h0P0eAhAUEQhDI1WSAAdgpALCAA4OC3wnKBxSkxuIn1e2M6uT42V5CHW3KJsvYvZa0H4gqhNz/OCg6uNghg7ffACRCOPJnDaF2T9qewnNTfsRwdBFEL6BtGEIQyShaIDmOl65WlQHGWtG3PYuv39f0/6fqYjwGNrlbDcxq5iBGvX/xbiM5oM0KYnghvIfSRCwiVChj7ibCed8FtQ3UJ8f/TGGy7H0G4CRIQBEEoo+SEN24ZMPR1Yb2yGMg9a38/Kg3Q62FhXesHdH/APWN0BrUG6CNKO93+NmG5OAs4vIJdFvs5iAWE0k252/1A70fZZXnIal3CMFILiq8ltyIaJCQgCIJQRi4gmltCFW95Euh4F7tcVQpkWnwFBjyvvJ8oWbSFWqPcry6I6yoshyQA7SwiYuubwMXd7LIhSOgTHCcsD5ylvE/OF+KvBcDJ39w3VmdYNRVY84SwftNU74yDaFSQgCAIQhmtQbo+QZQPgrvJVpYIT949Jyv7NQTLnPm8KSCCRIJApWKLhgHAjfNCu9jSIHaUDIhU3qc4G+cPD9Z+jK5w9Edhuct4xxxgCaKWUB4IgiCUET+JT1ghDVvkbppl+azXP8CmdA6KY6MwxDSXJVmyVS67LkjqBwz6t2AV8VcQBeLjFosDW2GgjtQDqUvEYbUE4UFIQBAEoYx/OHDvMkCjB9qNkm7jbpol2dK2sETr/bQdKV33poBQqYBBLwrr/uHWfcQ+EOLS3mJhIcbbAsIsq90hTjFOEB6EpjAIgrBNx7HW4gEQ/COO/iS0aXTSKQIAiOkMRLZml7nMk/JIDm9iUHCMFLdFtxeWbdXtkAuIS3trPy5n4BJHAYBax2bLJIg6gCwQBEE4D+8PwEjbxTfc589J/SgmrmETTHUZ7/HhOYxSZIW4LaE3cNfn0mgMOXJn0y+HAv/JrLtMkJyAUGmA/2ZT/geiziABQRCE89jKcth2FDBqPhDfw9rpMDhOGs7pCyhNS8jLc3cZZ38fSuW8c05LIz48CRe+aQgi8UDUKSQgCIJwnqBY5XaVCuhdj0IIlaYwnPXRUPKBuJ5WhwLCYoGw5aNBEB6C5CpBEM4j93Wor3CVNcXIw1drQilkMueMa+NxlMxjwOX97DJXw0PJEkIQHoQsEARBOE+AjZDG+ob8qb3dbc4X+QpuAiT2Y5fDkoBD33q2sJbZBCwZwQqH6buAEks1UVt5KgjCQ5AFgiAI15h10dsjqD3GUGE5MBaYsNx2tIUt1Brg4bXsX8vBbFu+B/83pTcEq8Ph74ViZvLqoAThYUhAEAThGsYQYNLv7LK4xkR9QqsXlt1RP4LbR85ptj6FJxDn3tj1IXA+lV0OjPHM5xGEDUhAEAThOs37s5aI+px74P4fWYfH2z+o/b5iOgFaI1ByHfj7k5r7u8L5bdL1M+vZVy4tN0HUESQgCIKoHcYQ583+vkSbYcCj24HodrXfl84opO7+awFQXVn7fcpZ/5JyO/lAEHUMCQiCIAh3cu9SNv13STZw7bB7921vfxTGSdQxJCAIgiDcid4faJXCLp9yc3nvtc/a+VwK4yTqFhIQBEEQ7oYTEAe/dd8+qyuBa0fYZaWy6d4u6kU0OkhAEARBuJu2lgJkZXnui8YovAKYKgCNARjwvPV2skAQdQwlkiIIgnA3XHZKxgyYKp3PbilmzQw28qLAklsiOA5WRcwAskAQdQ4JCIIgCHcjrtBZVVo7AXHwG+l6UBwrTOSQBYKoY2gKgyAIwt1odEJRrqoy1/ejFAaq8wMiWlu3G0hAEHWLxwTEm2++ib59+8Lf3x+hoaEOvYdhGLzyyiuIi4uDn58fUlJScOaMh4vSEARBeALOClFZ6vo+qkqs27o9AMR1AcZ9zZZO59D6uf45BOECHhMQlZWVuPfeezF9+nSH3/P222/jww8/xOLFi7Fnzx4EBARg+PDhKC8v99QwCYIgPAMnIC7tAf54HijJdX4fcvFxx0dA53vY5Q5jgJumAD0msk6VajIoE3WLx3wg5syZAwBYunSpQ/0ZhsH777+P//73vxgzZgwA4Ouvv0ZMTAxWr16NCRMmeGqoBEEQ7kdnsQiseZx9Lc0F7lni3D4qZRaI2M7SdZUKuGOha+MjiFriM5I1PT0dmZmZSElJ4dtCQkKQnJyM3bt323xfRUUFCgsLJX8EQRBeR+xICbiWlVI+hRHV1vXxEISb8RkBkZmZCQCIiZFWlIuJieG3KTF37lyEhITwfwkJCR4dJ0EQhEPo5D4JTtQLSVsHbJwNVBQJbUNeUdgnQXgPpwTErFmzoFKp7P6dOnXKU2NV5KWXXkJBQQH/d+nSpTr9fIIgCEXkN3tnCo6tGA/sfB/Y+zm7HtcV6G8njTVBeAGnfCCeffZZTJ482W6fFi1auDSQ2NhYAEBWVhbi4uL49qysLHTr1s3m+wwGAwyGWsRYEwRBeAJjiKzBhYql6ZbS3ZTjgfBBnBIQUVFRiIqK8shAmjdvjtjYWGzevJkXDIWFhdizZ49TkRwEQRA+gV+YdN2VkuflBeyrMbTWwyEId+MxH4iLFy/i0KFDuHjxIkwmEw4dOoRDhw6huLiY79OuXTv88ssvAACVSoWZM2fijTfewK+//oqjR49i4sSJiI+Px9ixYz01TIIgCM/gHyFrcFBAVFdYtzXtWevhEIS78VgY5yuvvIJly5bx6927dwcAbN26FYMGDQIApKWloaCggO/zwgsvoKSkBNOmTUN+fj769euHdevWwWg0emqYBEEQnsE/XLqulH5aCbHjJEeTXrUfD0G4GRXDuKtUnG9QWFiIkJAQFBQUIDg42NvDIQiisbLnU+DPF6RtT58AQprYf9+N88CH3Z1/H0G4AWfuoT4TxkkQBNGgkCd9AoAz62t+n5IFIijOuo0gvAxV4yQIgvAEzfoAD/wExHYB3m3Dtl09VPP7yi3J8AJjgKBYICGZ0lQTPgkJCIIgCE+gUgGth7LLdywEfv0/oDir5veV57OvoYnAlI0eGx5B1BaStQRBEJ7GEMS+Kk1PyCm9wb7KnTAJwscgAUEQBOFpeAHhQK2eMouA8CMBQfg2JCAIgiA8jcHize6MBUKeiIogfAwSEARBEJ6GS0XtiIAoy2Nf/UlAEL4NCQiCIAhP44wPBCcgyAJB+DgkIAiCIDwNJyBMlcqpqs9vAz7uC1zaC1SWsG36oLobH0G4AAkIgiAIT2MIAjR6dvnKAevt34wFso8DX48RBISBKnASvg0JCIIgCE+j1gDtRrPL57dab+fqZFSVApWWgoP6gLoZG0G4CAkIgiCIuiCiFftammu/Hy8gyAJB+DYkIAiCIOoCLq8DF6bJwU1ZyNfJAkH4OCQgCIIg6gIus2SZTEB8PkS6zlkoSEAQPg4JCIIgiLqAs0DkZUjbr59U7k9TGISPQwKCIAiiLuAsEHnpwLXD7DLD2O5PFgjCxyEBQRAEURfEdBSW937Ovsr9HzhCE4WwT4LwUUhAEARB1AU6P2DcN+zy9VPsqzyks0lPwBgKjP2ELQdOED6M1tsDIAiCaDQExrCvJdfZV84SwfGvDQAYQKOr02ERhCuQgCAIgqgrAiLZ17wLrDOlfApDQ5dkov5A31aCIIi6IiBKWP6gCxAY672xEEQtIR8IgiCIusIgK5BVnMm+BsYCd31u3Z8gfBiyQBAEQdQVSo6RfmHAs6fIaZKod5AFgiAIwpvEdSXxQNRLSEAQBEF4k7iu3h4BQbgECQiCIAhvEtfN2yMgCJcgAUEQBFGXJPaTrif1U+5HED4OOVESBEHUJfcsAfZ/BWiNQNuRQGC0t0dEEC5BAoIgCKIuCYoBBs3y9igIotbQFAZBEARBEE5DAoIgCIIgCKchAUEQBEEQhNOQgCAIgiAIwmlIQBAEQRAE4TQkIAiCIAiCcBoSEARBEARBOA0JCIIgCIIgnIYEBEEQBEEQTtPgMlEyDAMAKCws9PJICIIgCKJ+wd07uXupPRqcgCgqKgIAJCQkeHkkBEEQBFE/KSoqQkhIiN0+KsYRmVGPMJvNuHr1KoKCgqBSqdyyz8LCQiQkJODSpUsIDg52yz69DR1T/YCOqX7Q0I6poR0PQMfkKAzDoKioCPHx8VCr7Xs5NDgLhFqtRtOmTT2y7+Dg4AbzxeOgY6of0DHVDxraMTW04wHomByhJssDBzlREgRBEAThNCQgCIIgCIJwGhIQDmAwGDB79mwYDAZvD8Vt0DHVD+iY6gcN7Zga2vEAdEyeoME5URIEQRAE4XnIAkEQBEEQhNOQgCAIgiAIwmlIQBAEQRAE4TQkIAiCIAiCcJpGKyAWLVqEpKQkGI1GJCcnY+/evXb7//jjj2jXrh2MRiM6d+6MP/74Q7KdYRi88soriIuLg5+fH1JSUnDmzBlPHoIVzhzT559/jv79+yMsLAxhYWFISUmx6j958mSoVCrJ34gRIzx9GDzOHM/SpUutxmo0GiV96ts5GjRokNUxqVQqjB49mu/j7XO0fft23H777YiPj4dKpcLq1atrfE9qaip69OgBg8GAVq1aYenSpVZ9nP19uhNnj2nVqlUYOnQooqKiEBwcjD59+mD9+vWSPq+++qrVeWrXrp0Hj0KKs8eUmpqq+N3LzMyU9KtP50npt6JSqdCxY0e+jzfP09y5c3HTTTchKCgI0dHRGDt2LNLS0mp8nzfvTY1SQPzwww945plnMHv2bBw4cABdu3bF8OHDkZ2drdh/165duO+++/DII4/g4MGDGDt2LMaOHYtjx47xfd5++218+OGHWLx4Mfbs2YOAgAAMHz4c5eXlPnlMqampuO+++7B161bs3r0bCQkJGDZsGK5cuSLpN2LECFy7do3/W7FiRV0cjtPHA7DZ2MRjzcjIkGyvb+do1apVkuM5duwYNBoN7r33Xkk/b50jACgpKUHXrl2xaNEih/qnp6dj9OjRGDx4MA4dOoSZM2diypQpkhuuK+fenTh7TNu3b8fQoUPxxx9/YP/+/Rg8eDBuv/12HDx4UNKvY8eOkvP0119/eWL4ijh7TBxpaWmSMUdHR/Pb6tt5+uCDDyTHcunSJYSHh1v9nrx1nrZt24YnnngCf//9NzZu3IiqqioMGzYMJSUlNt/j9XsT0wjp3bs388QTT/DrJpOJiY+PZ+bOnavYf9y4cczo0aMlbcnJycyjjz7KMAzDmM1mJjY2lnnnnXf47fn5+YzBYGBWrFjhgSOwxtljklNdXc0EBQUxy5Yt49smTZrEjBkzxt1DdQhnj+err75iQkJCbO6vIZyjBQsWMEFBQUxxcTHf5s1zJAcA88svv9jt88ILLzAdO3aUtI0fP54ZPnw4v17b/5M7ceSYlOjQoQMzZ84cfn327NlM165d3TewWuDIMW3dupUBwOTl5dnsU9/P0y+//MKoVCrmwoULfJsvnafs7GwGALNt2zabfbx9b2p0FojKykrs378fKSkpfJtarUZKSgp2796t+J7du3dL+gPA8OHD+f7p6enIzMyU9AkJCUFycrLNfboTV45JTmlpKaqqqhAeHi5pT01NRXR0NNq2bYvp06cjNzfXrWNXwtXjKS4uRmJiIhISEjBmzBgcP36c39YQztGXX36JCRMmICAgQNLujXPkKjX9ltzxf/I2ZrMZRUVFVr+lM2fOID4+Hi1atMADDzyAixcvemmEjtOtWzfExcVh6NCh2LlzJ9/eEM7Tl19+iZSUFCQmJkrafeU8FRQUAIDV90iMt+9NjU5A5OTkwGQyISYmRtIeExNjNb/HkZmZabc/9+rMPt2JK8ck58UXX0R8fLzkizZixAh8/fXX2Lx5M9566y1s27YNI0eOhMlkcuv45bhyPG3btsWSJUuwZs0afPvttzCbzejbty8uX74MoP6fo7179+LYsWOYMmWKpN1b58hVbP2WCgsLUVZW5pbvsreZP38+iouLMW7cOL4tOTkZS5cuxbp16/DJJ58gPT0d/fv3R1FRkRdHapu4uDgsXrwYP//8M37++WckJCRg0KBBOHDgAAD3XHO8ydWrV/Hnn39a/Z585TyZzWbMnDkTt9xyCzp16mSzn7fvTQ2uGifhPPPmzcP333+P1NRUiePhhAkT+OXOnTujS5cuaNmyJVJTUzFkyBBvDNUmffr0QZ8+ffj1vn37on379vj000/x+uuve3Fk7uHLL79E586d0bt3b0l7fTpHjYHvvvsOc+bMwZo1ayT+AiNHjuSXu3TpguTkZCQmJmLlypV45JFHvDFUu7Rt2xZt27bl1/v27Ytz585hwYIF+Oabb7w4MvewbNkyhIaGYuzYsZJ2XzlPTzzxBI4dO1anfjKu0OgsEJGRkdBoNMjKypK0Z2VlITY2VvE9sbGxdvtzr87s0524ckwc8+fPx7x587BhwwZ06dLFbt8WLVogMjISZ8+erfWY7VGb4+HQ6XTo3r07P9b6fI5KSkrw/fffO3QBq6tz5Cq2fkvBwcHw8/Nzy7n3Ft9//z2mTJmClStXWpmV5YSGhqJNmzY+e56U6N27Nz/e+nyeGIbBkiVL8NBDD0Gv19vt643zNGPGDPz+++/YunUrmjZtarevt+9NjU5A6PV69OzZE5s3b+bbzGYzNm/eLHmCFdOnTx9JfwDYuHEj37958+aIjY2V9CksLMSePXts7tOduHJMAOud+/rrr2PdunXo1atXjZ9z+fJl5ObmIi4uzi3jtoWrxyPGZDLh6NGj/Fjr6zkC2DCtiooKPPjggzV+Tl2dI1ep6bfkjnPvDVasWIGHH34YK1askITZ2qK4uBjnzp3z2fOkxKFDh/jx1tfzBLDRDmfPnnVIkNfleWIYBjNmzMAvv/yCLVu2oHnz5jW+x+v3plq7YdZDvv/+e8ZgMDBLly5lTpw4wUybNo0JDQ1lMjMzGYZhmIceeoiZNWsW33/nzp2MVqtl5s+fz5w8eZKZPXs2o9PpmKNHj/J95s2bx4SGhjJr1qxhjhw5wowZM4Zp3rw5U1ZW5pPHNG/ePEav1zM//fQTc+3aNf6vqKiIYRiGKSoqYp577jlm9+7dTHp6OrNp0yamR48eTOvWrZny8nKfO545c+Yw69evZ86dO8fs37+fmTBhAmM0Gpnjx49Ljrk+nSOOfv36MePHj7dq9/Y54sZw8OBB5uDBgwwA5r333mMOHjzIZGRkMAzDMLNmzWIeeughvv/58+cZf39/5vnnn2dOnjzJLFq0iNFoNMy6dev4PjX9n3ztmJYvX85otVpm0aJFkt9Sfn4+3+fZZ59lUlNTmfT0dGbnzp1MSkoKExkZyWRnZ/vkMS1YsIBZvXo1c+bMGebo0aPMU089xajVambTpk18n/p2njgefPBBJjk5WXGf3jxP06dPZ0JCQpjU1FTJ96i0tJTv42v3pkYpIBiGYRYuXMg0a9aM0ev1TO/evZm///6b3zZw4EBm0qRJkv4rV65k2rRpw+j1eqZjx47M2rVrJdvNZjPz8ssvMzExMYzBYGCGDBnCpKWl1cWh8DhzTImJiQwAq7/Zs2czDMMwpaWlzLBhw5ioqChGp9MxiYmJzNSpU+vs4uDs8cycOZPvGxMTw4waNYo5cOCAZH/17RwxDMOcOnWKAcBs2LDBal++cI64cD/5H3cckyZNYgYOHGj1nm7dujF6vZ5p0aIF89VXX1nt197/ydM4e0wDBw60259h2FDVuLg4Rq/XM02aNGHGjx/PnD171meP6a233mJatmzJGI1GJjw8nBk0aBCzZcsWq/3Wp/PEMGwIo5+fH/PZZ58p7tOb50npWABIfh++dm+ict4EQRAEQThNo/OBIAiCIAii9pCAIAiCIAjCaUhAEARBEAThNCQgCIIgCIJwGhIQBEEQBEE4DQkIgiAIgiCchgQEQRAEQRBOQwKCIAiCIAinIQFBEARBEITTkIAgCIIgCMJpSEAQBEEQBOE0JCAIgiAIgnCa/wfsAz8qtrufRAAAAABJRU5ErkJggg==", + "text/plain": [ + "
" + ] + }, + "metadata": {}, + "output_type": "display_data" + } + ], + "source": [ + "sol = diffrax.diffeqsolve(\n", + " terms, diffrax.SPaRK(), t0, t1, dt0, y0, args, saveat=diffrax.SaveAt(steps=True)\n", + ")\n", + "\n", + "# Plotting the solution on ax1 and the BM on ax2\n", + "fig, (ax1, ax2) = plt.subplots(2, 1, figsize=(6, 8))\n", + "ax1.plot(sol.ts, sol.ys[:, 0], label=\"y_1\")\n", + "ax1.plot(sol.ts, sol.ys[:, 1], label=\"y_2\")\n", + "ax1.plot(sol.ts, sol.ys[:, 2], label=\"y_3\")\n", + "ax1.set_title(\"SDE solution\")\n", + "ax1.legend()\n", + "\n", + "bm_vals = jax.vmap(lambda t: bm.evaluate(t0, t))(jnp.clip(sol.ts, t0, t1))\n", + "ax2.plot(sol.ts, bm_vals[:, 0], label=\"BM_1\")\n", + "ax2.plot(sol.ts, bm_vals[:, 1], label=\"BM_2\")\n", + "ax2.set_title(\"Brownian motion\")\n", + "\n", + "plt.show()" + ] + }, + { + "cell_type": "markdown", + "id": "fd3251c814306cd", + "metadata": {}, + "source": [ + "## Using adaptive time-stepping via the PID-controller\n", + "\n", + "In order to use adaptive time stepping, the solver must produce an estimate of its error on each step. This is then used by the PID controller to adjust the step size.\n", + "To perform this error estimation the `SPaRK` solver uses an embedded method. For solvers like `GeneralShARK`, which do not have an embedded method, we'd instead need to use `HalfSolver(GeneralShARK())` as the solver in order to estimate the error." + ] + }, + { + "cell_type": "code", + "execution_count": 5, + "id": "42ca5c5520079b5f", + "metadata": { + "ExecuteTime": { + "end_time": "2025-10-09T01:03:43.605326Z", + "start_time": "2025-10-09T01:03:35.705678Z" + } + }, + "outputs": [ + { + "name": "stdout", + "output_type": "stream", + "text": [ + "Accepted steps: 2968, Rejected steps: 1637, total steps: 4605\n" + ] + }, + { + "data": { + "image/png": "iVBORw0KGgoAAAANSUhEUgAAAgsAAAKqCAYAAAC0Hx2FAAAAOXRFWHRTb2Z0d2FyZQBNYXRwbG90bGliIHZlcnNpb24zLjkuMiwgaHR0cHM6Ly9tYXRwbG90bGliLm9yZy8hTgPZAAAACXBIWXMAAA9hAAAPYQGoP6dpAADIGUlEQVR4nOydd1xT1/vHP0nYe4MIghMUERUX7o2jVq21rVbrbmvV1tpph6OtWkcdbR2tP+v4Wne17l1xb8W9BUUUUNkrCcn9/XHJuMnNgoQEfN6vV17ee8655z5JMOe5z3mGgGEYBgRBEARBEDoQWlsAgiAIgiBsG1IWCIIgCILQCykLBEEQBEHohZQFgiAIgiD0QsoCQRAEQRB6IWWBIAiCIAi9kLJAEARBEIReSFkgCIIgCEIvpCwQBEEQBKEXUhYIgrAa06ZNg0AgMPu84eHhGD58uNnnJYhXFVIWCMJGuXbtGt58802EhYXByckJ1atXR7du3fDbb79xxoWHh0MgEEAgEEAoFMLLywvR0dF4//33cfbsWd65FeP5Xh9++GFFvL1yc+rUKUybNg3Z2dnWFoUgqjx21haAIAhtTp06hU6dOqFGjRoYM2YMgoKCkJKSgjNnzmDRokWYMGECZ3zjxo3x2WefAQDy8vJw69YtbN68GcuXL8enn36K+fPna92jW7dueO+997Ta69WrZ5k3ZWZOnTqF6dOnY/jw4fDy8uL03blzB0IhPQsRhLkgZYEgbJAZM2bA09MT58+f11oIMzIytMZXr14dQ4YM4bTNnj0bgwcPxoIFC1C3bl2MHTuW01+vXj2ta6oKjo6O1haBIKoUpHoThA3y4MEDREVFaSkKABAQEGDUHM7Ozvjf//4HHx8fzJgxA+YqMJuXl4eJEyciPDwcjo6OCAgIQLdu3XDp0iXOuM2bNyM2NhbOzs7w8/PDkCFDkJqaqnfu5ORkCAQCrFq1SqtPIBBg2rRpAFhfhy+++AIAULNmTeUWSnJyMgB+n4WHDx9i4MCB8PHxgYuLC1q1aoXdu3dzxiQkJEAgEGDTpk2YMWMGQkJC4OTkhC5duuD+/fvGf0gEUcUgywJB2CBhYWE4ffo0rl+/joYNG5Z5Hjc3N/Tv3x8rVqzAzZs3ERUVpewrLi7GixcvtK7x8PCAg4ODzjk//PBDbNmyBePHj0eDBg3w8uVLnDhxArdu3ULTpk0BAKtWrcKIESPQvHlzzJo1C+np6Vi0aBFOnjyJy5cv8ypBpvDGG2/g7t27WL9+PRYsWAA/Pz8AgL+/P+/49PR0tG7dGoWFhfj444/h6+uL1atX4/XXX8eWLVvQv39/zviff/4ZQqEQn3/+OXJycjBnzhy8++67On1ACKLKwxAEYXMcOHCAEYlEjEgkYuLi4pgvv/yS2b9/PyORSLTGhoWFMb1799Y514IFCxgAzPbt25VtAHS+1q9fr1c2T09PZty4cTr7JRIJExAQwDRs2JApKipStu/atYsBwEyZMkXZNnXqVEb9ZygpKYkBwKxcuVJrXgDM1KlTledz585lADBJSUlaY8PCwphhw4YpzydOnMgAYI4fP65sy8vLY2rWrMmEh4czMpmMYRiGOXLkCAOAqV+/PiMWi5VjFy1axABgrl27pvN9E0RVhrYhCMIG6datG06fPo3XX38dV65cwZw5cxAfH4/q1atjx44dJs3l5uYGgN0+UKdv3744ePCg1qtTp0565/Py8sLZs2fx9OlT3v4LFy4gIyMDH330EZycnJTtvXv3RmRkpJbpvyLYs2cPWrRogbZt2yrb3Nzc8P777yM5ORk3b97kjB8xYgTHutKuXTsA7FYGQbyK0DYEQdgozZs3x9atWyGRSHDlyhVs27YNCxYswJtvvonExEQ0aNDAqHny8/MBAO7u7pz2kJAQdO3a1WS55syZg2HDhiE0NBSxsbHo1asX3nvvPdSqVQsA8OjRIwBARESE1rWRkZE4ceKEyfcsL48ePULLli212uvXr6/sV9/uqVGjBmect7c3ACArK8uCUhKE7UKWBYKwcRwcHNC8eXPMnDkTS5cuhVQqxebNm42+/vr16wCAOnXqmEWet956Cw8fPsRvv/2G4OBgzJ07F1FRUdi7d2+559aVoEkmk5V7blMQiUS87YyZnEQJorJBygJBVCKaNWsGAHj27JlR4/Pz87Ft2zaEhoYqn6LNQbVq1fDRRx/h33//RVJSEnx9fTFjxgwArHMmwOY60OTOnTvKfj4UT/CaiZYU1gp1TMn8GBYWxivP7du3OTITBMEPKQsEYYMcOXKE9yl2z549APhN/JoUFRVh6NChyMzMxLfffmuWtMoymQw5OTmctoCAAAQHB0MsFgNgFZqAgAAsW7ZM2QYAe/fuxa1bt9C7d2+d83t4eMDPzw/Hjh3jtC9ZskRrrKurKwBtxYKPXr164dy5czh9+rSyraCgAH/++SfCw8ON3tIhiFcV8lkgCBtkwoQJKCwsRP/+/REZGQmJRIJTp05h48aNCA8Px4gRIzjjU1NTsXbtWgCsNeHmzZvYvHkz0tLS8Nlnn+GDDz7Qusfdu3eV16gTGBiIbt268cqVl5eHkJAQvPnmm4iJiYGbmxsOHTqE8+fP45dffgEA2NvbY/bs2RgxYgQ6dOiAQYMGKUMnw8PD8emnn+p976NHj8bPP/+M0aNHo1mzZjh27Bju3r2rNS42NhYA8O233+Kdd96Bvb09+vTpo1Qi1Pn666+xfv169OzZEx9//DF8fHywevVqJCUl4Z9//qFsjwRhCGuHYxAEoc3evXuZkSNHMpGRkYybmxvj4ODA1KlTh5kwYQKTnp7OGRsWFqYMexQIBIyHhwcTFRXFjBkzhjl79izv/NATOtmhQwedconFYuaLL75gYmJiGHd3d8bV1ZWJiYlhlixZojV248aNTJMmTRhHR0fGx8eHeffdd5knT55wxmiGTjIMwxQWFjKjRo1iPD09GXd3d+att95iMjIytEInGYZhfvzxR6Z69eqMUCjkhFFqhk4yDMM8ePCAefPNNxkvLy/GycmJadGiBbNr1y7OGEXo5ObNmznt+kI6CeJVQMAw5LFDEARBEIRuyPZGEARBEIReSFkgCIIgCEIvpCwQBEEQBKEXUhYIgiAIgtCLRZWFadOmKUvHKl6RkZHK/uLiYowbNw6+vr5wc3PDgAEDkJ6ebkmRCIIgCIIwEYtbFqKiovDs2TPlSz0v/KeffoqdO3di8+bNOHr0KJ4+fYo33njD0iIRBEEQBGECFk/KZGdnh6CgIK32nJwcrFixAuvWrUPnzp0BACtXrkT9+vVx5swZtGrVyqj55XI5nj59Cnd3d7NkqCMIgiCIVwWGYZCXl4fg4GC9ycksrizcu3cPwcHBcHJyQlxcHGbNmoUaNWrg4sWLkEqlnKp3kZGRqFGjBk6fPm20svD06VOEhoZaSnyCIAiCqPKkpKQgJCREZ79FlYWWLVti1apViIiIwLNnzzB9+nS0a9cO169fR1paGhwcHODl5cW5JjAwEGlpaTrnFIvFnHzzipxSKSkp8PDwsMj7IAiCIIiqSG5uLkJDQ7VK2GtiUWWhZ8+eyuNGjRqhZcuWCAsLw6ZNm+Ds7FymOWfNmoXp06drtXt4eJCyQBAEQRBlwNA2foWGTnp5eaFevXq4f/8+goKCIJFItCrGpaen8/o4KJg8eTJycnKUr5SUFAtLTRAEQRCvNhWqLOTn5+PBgweoVq0aYmNjYW9vj8OHDyv779y5g8ePHyMuLk7nHI6OjkorAlkTCIIgCMLyWHQb4vPPP0efPn0QFhaGp0+fYurUqRCJRBg0aBA8PT0xatQoTJo0CT4+PvDw8MCECRMQFxdntHMjQRAEQRCWx6LKwpMnTzBo0CC8fPkS/v7+aNu2Lc6cOQN/f38AwIIFCyAUCjFgwACIxWLEx8djyZIlFpFFJpNBKpVaZO7KjL29PUQikbXFIAiCIGyYSl+iOjc3F56ensjJyeHdkmAYBmlpaVq+EYQKLy8vBAUFUZ4KgiCIVwxDa6gCi+dZsDYKRSEgIAAuLi60IKrBMAwKCwuRkZEBAKhWrZqVJSIIgiBskSqtLMhkMqWi4Ovra21xbBJFCGtGRgYCAgJoS4IgCILQokpXnVT4KLi4uFhZEttG8fmQTwdBEATBR5VWFhTQ1oN+6PMhCIKwLWzNnfCVUBYIgiAIorKQVZyFrpu7YtbZWdYWRQkpCwRBEARhQ0xKmISMogysu73O2qIoIWWhivHs2TMMHjwY9erVg1AoxMSJE60tEkEQBGEEDMNg8vHJuJB+wdqiaEHKQhVDLBbD398f3333HWJiYqwtDkEQBGEkj/MeY9fDXdYWgxdSFmyQNWvWwNfXl1OKGwD69euHoUOH6r02PDwcixYtwnvvvQdPT09LikkQBEGYETkj12p7UfTCCpJo80opCwzDoFBSYpWXKZ6tAwcOhEwmw44dO5RtGRkZ2L17N0aOHGmJj4YgCIKwMgy014nLGZetIIk2VTopkyZFUhkaTNlvlXvf/CEeLg7GfdzOzs4YPHgwVq5ciYEDBwIA1q5dixo1aqBjx44WlJIgCIKwFnwPlTde3EC3sG5WkIbLK2VZqEyMGTMGBw4cQGpqKgBg1apVGD58OOVEIAiCqKLIGJlW24rrK6wgiTavlGXB2V6Emz/EW+3eptCkSRPExMRgzZo16N69O27cuIHdu3dbSDqCIAjC2thaIiZ1XillQSAQGL0VYAuMHj0aCxcuRGpqKrp27YrQ0FBri0QQBEFYiKScJGuLoBPahrBhBg8ejCdPnmD58uUmOTYmJiYiMTER+fn5eP78ORITE3Hz5k0LSkoQBEGUly+OfWFtEXRSeR6zX0E8PT0xYMAA7N69G/369TP6uiZNmiiPL168iHXr1iEsLAzJycnmF5IgCIKo8pBlwcZJTU3Fu+++C0dHR6OvYRhG60WKAkEQhO2RWZyJJmuaYNu9bTrH5IhzKlAifkhZsFGysrKwbds2JCQkYNy4cdYWhyAIgrAAHTZ2QAlTgimnpugcYwu5FkhZsFGaNGmC4cOHY/bs2YiIiFC2R0VFwc3Njff1999/W1FigiAIwlzU9qytPL6VecuKkrCQz4KNomvbYM+ePZBKpbx9gYGBFpSIIAiCqCgmNJmAiQkTAQBLEpdgbMxYq8pDykIlIywszNoiEARBEGaikV8jXH1xVbvDxvLv0TYEQRAEQViJOt51eNsFNqYtkLJAEARBEDZE39p9OecBzgFWkkQFKQsEQRAEYSX4Ujx/2+pbzrmHo0dFiaMTUhYIgiAI2+DlA2BBQ+Dsn9aWpMLgK0vtbOfM2YawhZoRFaYs/PzzzxAIBJg4caKyrbi4GOPGjYOvry/c3NwwYMAApKenV5RIBEEQhC2x72sgJwXYa7tpj81NgbTA4JgSpqQCJNFPhSgL58+fxx9//IFGjRpx2j/99FPs3LkTmzdvxtGjR/H06VO88cYbFSESQRAEYWuUiK0tQYWRWZyJWy9v4eCjgwbHlshfAWUhPz8f7777LpYvXw5vb29le05ODlasWIH58+ejc+fOiI2NxcqVK3Hq1CmcOXPG0mJVWbZu3Ypu3brB398fHh4eiIuLw/79+60tFkEQhOnkPgVk/HllKjN/3/obHTZ2wFu73tLq+7rF11ptUjn7GRRKCyGTyywuHx8WVxbGjRuH3r17o2vXrpz2ixcvQiqVctojIyNRo0YNnD592tJiVVmOHTuGbt26Yc+ePbh48SI6deqEPn364PJl66cLJQiC0EvSUdXxpveA+fWBFd3NMzfDAKteA3Z+Yp75ysHP537W2edi5wIAiPBRZe4tkZcgozADLde1xMj9xlcgNicWVRY2bNiAS5cuYdasWVp9aWlpcHBwgJeXF6c9MDAQaWlpOucUi8XIzc3lvKoaa9asga+vL8RirkmuX79+GDp0qN5rFy5ciC+//BLNmzdH3bp1MXPmTNStWxc7d+60pMgEQRDm5eZ29t+nl8wz3529QPJx4OIqm7ZWCASsY2OwWzBmt5sNgFUW9iezFuJLGWb6PEzEYspCSkoKPvnkE/z9999wcnIy27yzZs2Cp6en8hUaGmr8xQwDSAqs8zLBm3XgwIGQyWTYsWOHsi0jIwO7d+/GyJGmaZVyuRx5eXnw8fEx6TqCIIgqwY6Pgb96AvfVfAMSbbeOjlCgWpaj/aIBsMpCgIt1cy1YLN3zxYsXkZGRgaZNmyrbZDIZjh07ht9//x379++HRCJBdnY2x7qQnp6OoKAgnfNOnjwZkyZNUp7n5uYarzBIC4GZwSa/F7PwzVPAwdWooc7Ozhg8eDBWrlyJgQMHAgDWrl2LGjVqoGPHjibddt68ecjPz8dbb2nvjREEQVR5Lq1m/3X1VbVd3wrEDq9wUeSMHBKZBEKBEHJGzjumRVAL5bGdkF2ipXIp3OzdAACRPpGWF5QHiykLXbp0wbVr1zhtI0aMQGRkJL766iuEhobC3t4ehw8fxoABAwAAd+7cwePHjxEXF6dzXkdHRzg6OlpKbJthzJgxaN68OVJTU1G9enWsWrUKw4cPV5qojGHdunWYPn06tm/fjoAA62cAIwiC0IlcY/Hs/wew7YPyzTnNU3V8S20rVt03ogKQyWX489qfWJK4BCKBSKeiAABBrqqHZYWyUCIvUTo52gvtLSusDiymLLi7u6Nhw4acNldXV/j6+irbR40ahUmTJsHHxwceHh6YMGEC4uLi0KpVK8sIZe/CPuFbA3sXk4Y3adIEMTExWLNmDbp3744bN25g9+7dRl+/YcMGjB49Gps3b9ZyLiUIgrA5ZGo+Wh8nAqLSRbGsi+ORmdzzoEZAGk/BpgpgT9IeLElcAgCQMcZHMziIHACwiZuKSooAVEFlwRgWLFgAoVCIAQMGQCwWIz4+HkuWLLHcDQUCo7cCbIHRo0dj4cKFSE1NRdeuXY3eblm/fj1GjhyJDRs2oHfv3haWkiAIwgyUFKuOPUMAcR57LJeyDokiExfJo7O555qKAsOwa0IFkJKXYtQ4L0cvzrmjSGVFz5Own4e9qZ+DmahQZSEhIYFz7uTkhMWLF2Px4sUVKUalYfDgwfj888+xfPlyrFmzxqhr1q1bh2HDhmHRokVo2bKlMrLE2dkZnp6eBq4mCIKwEoqETAIhILTjPthJCgBnL/Pe78VdwD/C8DgzYOz2sWLbQQGvsmAlywLVhrBhPD09MWDAALi5uaFfv35GXfPnn3+ipKQE48aNQ7Vq1ZSvTz6xfmwxQRCEThTKgp0T+8QvcgAEIrZNWmj++y1uAWQ9Mv+8GjAMg50PjAtd11QWBAIBnO2cAVhfWbDqNgRhmNTUVLz77rtGO3VqWm8IgiAqBUplofS3TrFtLM4FJCYqCwlqWxBhbYBHJ/nHLWoETE4FHN1Ml9dIbry8YfQ2RC3PWlptznbOKCopwsOchwBUfgwVDVkWbJSsrCxs27YNCQkJGDdunLXFIQiCYCmRmJQ3xvh5S30W1Ezvyq2I32OBKxuNnytBzbmx4QD9Yy+sMH7eMnAv657BMXPbz0XPmj3xY5sftfpc7dnP4EjKEQC0DUFo0KRJEwwfPhyzZ89GRIRqXy0qKgpubm68r7//tt1EIwRBVAGKsoGf/IGlrdlzY5WG5JPAiQWATE9BpKulykC+WgZfB7Un/m3vG74Pw3DDJQHA1Z973kPD8bE4R+VMaSXiguMwp/0c3sRLivTPCmgbguCQnJzM275nzx5IpfypSgMDAy0oEUEQrzxza7P/ZtwEtowCXtwBRv8H2GmYxg//wC70bT9lz1f1Yv89NI39d1qO9twXV2u36YteO/Ubu9B3+lYV1bC6j+E5Gr3FKgdHfmLPj//Cvj46CwSYP+GRvpwKClz0hNYrLAsKrLUNQcpCJSMsLMzaIhAE8aqiXir5+hb23/2Tgd6/qNoLXrCLLwAcns4/T8Zt7YW5Xjw7Z8wgVduzRO6Yh0eBWh3YapQHvmPbanUEwtuyx8nHte+lqSw4eQEdvlApCwq2jAA+Mn8RQzkMKwv6rAXO9s5Gj7UktA1BEARBlJ3z/8c9f3LB8DX7NMowv3ygUj4UERB8rHmd3WpIOatqW1WaS0ZX6WYvjQcsYemyV68Htz3jpn6Zy0iJXM/WixFobkMoMjlWNKQsEARBEIYR5+vuKy6t/vvoNLD+bcNz+Wh4/R+dozq+/o/quNE72tdO9wI2D+e27foU+EGtWF6/papjj2qqY0cP1XHn77XnZhggL51NDa1L+TARkQ7lp0d4D952TRQ1IRRkFGaUW6ayQMoCQRAEYRj1DIua3NnL/rtSzwLoFgTU6sQe+9bh9oW3UR33VHNA7LsYGH8RmJqtX7YLf3HPg5vyj/OorjoOqK/dP90L+KUesHEIcPl/+u9pJGeeneFtn9BkAra+vhVnB5/l7VdwPu085zyrOMsscpkKKQsEQRCEYfQ56hkTqTBoPeBW6oStaUrfpaokDJ+aqmORHeBXx/S0zAGRwLBdrKLBQS16QygCBq7SPceNf027pw4OPjrI224vtEdd77p6nRsB4En+E8554vNEs8hlKqQsEARBEIZ5mKC/X7EVoQv3IHbxB4CDU4AX99njEwu1lQdjePcfIG68dvtXyey/NduxioY6mgpPVH92/MeXtecpp6+BPkY1HIVqbtUMDwTg4+RjeFAFQMoCQRAEYRhDORXUfQ34cA0A7NQ8+3+PBR4cAQ5N5Y4LbsJ//fsJ7L9v/Y8NvazbFYifAdTppjbmKODsrVsGAc+S5+zN+lD0/4PbboHcCxte24Bz757DxNiJRl/zVfOvOOd9a/c1s1TGQcpCFePEiRNo06YNfH194ezsjMjISCxYsMDaYhEEUdnRfErXgkeZUDgRugezVoVaHbj9/+vHPR+wAnB0558+uAmrJDR4nduefl11XJ6qwjHvAJ5qlX2fJbIOm+UgsziTcx7lG6Ws9WAsbaq34ZyLhHqiRSwI5VmoYri6umL8+PFo1KgRXF1dceLECXzwwQdwdXXF++8bsa9IEATBh6HogKNzueeDNgARPYGgRoBHMNtWr6f+OcqyEOY9Ux2rOzCq02wUm9aZLwJCnTafAHs+V52fXQaExZkuUykFkoIyX6vA09ETb9R9A1vvbQWgHR1RUZBlwQZZs2YNfH19IRaLOe39+vXD0KFD9V7bpEkTDBo0CFFRUQgPD8eQIUMQHx+P48d5kpUQBEHo4+hc1r8A4N/Dd/JSHXvV4PYpLAT1ugNBDdljkYHnU33poHXR6TvVsYMOZ8HevwCf3wPqv6Z/ruajgT6/qnI9JB0zXR41hELVEmso6kEf01tPx8y2MxFXLQ7vN7LOQ98rZVlgGAZFJUVWubeznbPRNc0HDhyIjz/+GDt27MDAgQMBABkZGdi9ezcOHDhg0n0vX76MU6dO4aeffjI8mCAIQoGsRJXlMHY49wlewZdJ7NP3/smAq5+qXWgPBDbknze8HX+mRQDwq2u6nB2+YB0VvUJ1jxEIADftugu842KHscrN8s5AUSYgKSjz9oas1Brjau9qMOrBEH1q90Gf2jzprCuIV0pZKCopQst1La1y77ODzxr9x+Ls7IzBgwdj5cqVSmVh7dq1qFGjBjp27GjUHCEhIXj+/DlKSkowbdo0jB49uqyiEwTxKqJe0OlXHU6HQqGqvHPBC1X71491P+V3mQqs6Krd/vZaILhxmUQ17E9hIt5q4ZsHvgNeK5vfVwnDWkp0JWaqTNA2hI0yZswYHDhwAKmpqQCAVatWYfjw4UZbJ44fP44LFy5g2bJlWLhwIdavX29JcQmCqGrs/lx/vyILo+Kpu7BUWRAIAXs9Tny6lIj61ntq1sJFLVxRM+GTCTwvfA4AyJUYCCutBLxSlgVnO+dy7RuV996m0KRJE8TExGDNmjXo3r07bty4gd27dxt9fc2arGYcHR2N9PR0TJs2DYMGDTJwFUEQRCl39/K3O3kCrT5ityYAwKHUN0FhWbBz0p9Eic+kP4o/cZFVCWsLPDqhPxTTALsfGv+bbeu8UsqCQCAo975RRTJ69GgsXLgQqamp6Nq1K0JD9ezJ6UEul2s5SxIEQZSJ4hygo1ohKMU2RHE2+y9fLgN11EsuN+jHZlE0NUNjRdBvMbAoBijKYiNByhCpUde7DD4YNgptQ9gwgwcPxpMnT7B8+XKMHDnSqGsWL16MnTt34t69e7h37x5WrFiBefPmYciQIRaWliCISg/DANM82ZexaD6ASfQUnAK4loWb/9qmogCwtSwU/Pdjmaa49vwaAKB7WHdzSGRVXinLQmXD09MTAwYMwO7du9GvXz+jrpHL5Zg8eTKSkpJgZ2eH2rVrY/bs2fjggw8sKyxBEJWftGumXyNyMG28Pn8GW8LeSXV8otTBses0oy+/9vwa9iazWzn2InszCmYdSFmwcVJTU/Huu+/C0dHRqPETJkzAhAkTLCwVQRBVEwMpnfkwdSG0VUsCHw0HqNJYn1gAxE0AXH2NulS94JOD0ESFygahbQgbJSsrC9u2bUNCQgLGjRtnbXEIgngVEJbh+VHzGnfjCiQBAPotM/1+FYlnCPd8bi2jL7VT+1zshZXfskDKgo3SpEkTDB8+HLNnz0ZERISyPSoqCm5ubryvv//+24oSEwRR6THknMiHpmWh31LD14w5wjo2NrbxCK0WH7DRHeo8PmPUpTPPzlQe0zaEAZYuXYqlS5ciOTkZALvQTZkyBT17svnBi4uL8dlnn2HDhg0Qi8WIj4/HkiVLEBgYaEmxKgWKz0yTPXv2QCrlL+dKnxtBEGUmYTaQMNPwOE00HRxrdzJ8TfWm7MvW8azOZqnMSgKWtmbb/opnC1rpQa5RCrsqbENYVFkICQnBzz//jLp164JhGKxevRp9+/bF5cuXERUVhU8//RS7d+/G5s2b4enpifHjx+ONN97AyZMnLSlWpSYsLMzaIhAEUdUQ55VNUQC4CYyqIg4uQGAUm75aUeGy4AU3vbUGUjn3gU4sq/yh6xbdhujTpw969eqFunXrol69epgxYwbc3Nxw5swZ5OTkYMWKFZg/fz46d+6M2NhYrFy5EqdOncKZM8aZeQiCIAgzcO5P/vYG/Yy7vntpDYlgHWmhqwKjD6uO59Zmw0sfHOEdKpVxlYViWbElJasQKsxnQSaTYcOGDSgoKEBcXBwuXrwIqVSKrl1VOcIjIyNRo0YNnD5dvhrimjBMGTx8XyHo8yGIV5wb/3LPQ5oDHxwD+v+hPbbTt9ptrT4Chv4LvLfDEtLZBvZO7Oeizv/68Q7VtCzYCSp/4KHFlYVr167Bzc0Njo6O+PDDD7Ft2zY0aNAAaWlpcHBwgJeXF2d8YGAg0tLS+CcDIBaLkZuby3npwt6edSopLCw0y3upqig+H8XnRRDEK4am38GT80C1GG6uAQB4bzvQ/gvt64Ui1lfBycNyMtoCIw8ANVobHKZZC+JICr8FojJhcXUnIiICiYmJyMnJwZYtWzBs2DAcPXq0zPPNmjUL06dPN2qsSCSCl5cXMjIyAAAuLi5GF2J6FWAYBoWFhcjIyICXlxdEospfGY0giDLgHqS7r/sM4ECpNcG3buXKk2BuhEJg5F4g9SJbwhpgS3lLC4Cdn7Dlt5uNxLIr3JDQ50XPrSCsebG4suDg4IA6ddjyobGxsTh//jwWLVqEt99+GxKJBNnZ2RzrQnp6OoKCdP/hTp48GZMmTVKe5+bm6q2ZoJhLoTAQ2nh5een9zAmCqOI4uuvu81Zzqi5LHoaqSDU134wf1ZI03diG60IZdj3cxRneyK9RBQlmOSr8m1cUNYqNjYW9vT0OHz6MAQMGAADu3LmDx48fIy4uTuf1jo6ORmczBNjiUdWqVUNAQIDOkMNXGXt7e7IoEMSrjlymu0+cpzomZYFFyL+Df8LZCWOv/6bVvrDTQgsLZHks+s1PnjwZPXv2RI0aNZCXl4d169YhISEB+/fvh6enJ0aNGoVJkybBx8cHHh4emDBhAuLi4tCqVSuzyyISiWhRJAiC4ENewj1//XfV8YP/VMc6FslXknafA8fncZpm+fKXs/Z38a8IiSyKRZWFjIwMvPfee3j27Bk8PT3RqFEj7N+/H926dQMALFiwAEKhEAMGDOAkZSIIgiAqEE1loelQ1bF6tJRdJSkCVRF0+R4IigYOTQVaTwDu7MNj2R2tYTXca1hBOPNjUWVhxYoVevudnJywePFiLF682JJiEARBEPrQVBbUqdURuL6FPdaMjjAzWcVZ+OH0D+hftz/ah7S36L3MQlQ/9gUA55YDbtpDmLIU57JBaAOKIAjiVeby38AtPfkRGg8GclKA8LYWF+WXC7/g0ONDOPT4EK4NU5XLzirOQnJuMhr7N7bdiDYnLwBVN0yfNqAIgiBeVfIzgO0f6R8jFAGdvgFqWv5JP61QlWMnR6yqvzBw50C8t/c9XEy/aHEZyoqs1Ye87c5VZOuGlAWCIIhXFfVIBwXVYipeDgVqFvsliSr/tfTCdADA4ceHNa+wGcS1Omi1OYucMLNtGWtu2Bi0DUEQBEGwdJkKtBprbSkAAE8LniJHnIPEjERlm4PIdqs33s5J0mrbH78G3j4RVpDG/JCyQBAE8apydz/33DscsLee2bywRLXnL4QQI/aPwL2se8o2J5FlHSzLwx/XtItx2Re+tIIkloG2IQiCIF5FnlwE9k/mtlk56dK1FyqnRpFQxFEUANu2LGQVZ2m1OazsDdzaxTO68kHKAkEQxKvIvzzbDTaeoXHhpYVIyU2x+H1K9IWS6qCopEirzQ5Q1dWo5JCyQBAE8SoiKdBuE9p+ltte23qBYSyXu+DPq3+i7Ya2uJ9136Tr+JQFdoG10VBPEyFlgSAI4lWEZ3HTWyPChjjw6IDF5v7t8m8okBZg3oV5hger0bNmT+XxrpSnOPg4lT2x1bwQJkLKAkEQxKsIn6ldnFvxcpSiaS1wtnNGhDd/JMHepL2Wl8fIzIuyUgVLJGCtMkNychFWUoIgmULxqhrKgm1vUBEEQRCWoThHu41va6KCKGG4ysuOB7qzSlZEvgVDWx3P8p+h+z/dAQA/tfkJ+5L3AQCcXfyBzGzVwCpiWSBlgSAIgmCJ6Gl4jIWQyqRWuzcfj3If6e1XKAoA8N3J75THzyO6A0/UojhyUs0umzWgbQiCIAiCxSPYareWym1LWXha8JQ3HNIQ+Zq+IHy+IZUQUhYIgiBeNTRN7K0+Aj69aR1ZSrE1ZQEA7mRpl5w2hKOdI1A33gLSWBdSFgiCIF411J0ba3cGus8APKtbTx6YntvAkuGTCvIkPLUzDOAkcgIYObfxhWlhmLYIKQsEQRCvGiXFquN31gFCyy4Fhx4dws/nflZGDvBhqs9CWRInGUJTASmLsuAgcgC6fM9tXDewPGLZBOTgSBAE8apxRK0SosjRYrdZnLgYPk4+mHmWvV+ASwBGNhzJO9bUbYjHeY9R26t2uWVU50neE845X6IlQL9V4+2ItwGRG7cx82G5ZbM2ZFkgCIJ41TijKv9sKatCUk4Sll1ZplQUAGDBxQU6xyuUBT9nP6PmFwrML/eE/yZwznUpCxK5ROcctb1qa6fN9o8st2zWhpQFgiAIwuy8LOKvuNjjnx687YqFmW+r4vy75xHmEYbuYd2VyoRYJjaTpFDe90HOA05bobSQdyzfvTuHdsaRt46wJyINZaFWJ7PIaE1IWSAIgiDMzuWMy7ztqfmpvFsOv1/+HQCQJc7CV82/4vQ52TlhR78d+KXjL3As3TYxp7KQI85B3Po4rfZ8aT7nPLM4E4N3D8bqG6u1xo5oOEJlFRHaczvVfUQqKaQsEARBEGanhkcNnX03XtzQaruXrUpk1KOmtvVBse2Qms8mOXpW8Ky8IirZ/XA375ZDtjibc/6/m//DtRfX8OfVP7XGcq4XaSgLzt7mENOqkLJAEATxqlJD+2naXOiLVlicuFirrXsYmxFxQN0BRvktfHH0i7ILp0GgSyBv++3M25xzfY6NDf0aqk40fRY0lYdKCEVDEARBvEqoe+b3nm+x20hkup0A+Z7iN9zZAACw01xoKwB7HYt5Uk4SZHIZRKWluz0dPbXGXBp6CeISMdwc1CIgBALAvz7w/BZ7ruezqCyQZYEgCOJVYo/aE7mLj9mnZxgGH//3MaacmqJzTPOg5pxzxdYCAGy8s9HsMhlCn2KjbgXxcPDQ6rcX2nMVBQXNR6mOS0hZIAiCICoT9w+pjvkWuXKSXpiOIylH9I7xdfLlnP966Vfl8YcxH5br/jniHCy/upyjgBhCn7Kw4vqKsgmiHtVh5sgNa2BRZWHWrFlo3rw53N3dERAQgH79+uHOHW6u7eLiYowbNw6+vr5wc3PDgAEDkJ6ebkmxCIIgCABwcDX7lLmSXINjZAw3PLJJQBPlcYugFgCA3zr/BgCY2XYmZ+zbEW/rnfu7k9/h18u/6gzR5ENf3gS5Wupmk7JGOqlZIUpIWdDL0aNHMW7cOJw5cwYHDx6EVCpF9+7dUVCgqpn+6aefYufOndi8eTOOHj2Kp0+f4o033rCkWARBEFWfghfAnb3cJ1yphq+AQGDytIceHcLV51dxJ/MOzjw7o9V/OZ0/ZFIdTWXBkSeLZMfQjrg27Br61O7DaY8PZ4s01fbkz96YkJJg8P6aZBdnGzWuhDFBWYhWS/FcBXwWLOpJsm/fPs75qlWrEBAQgIsXL6J9+/bIycnBihUrsG7dOnTu3BkAsHLlStSvXx9nzpxBq1atLCkeQRBE1WVu6WJasz0wbCd7LFardcDjrGeIfcn7tKIQtvfbjlqetQAAW+9txU9nf+K99uMmH+NhzkPsergLa26sQZ9afeDv4g8AvKGIurAvzWHAl2ehrMWlfrn4C+d8S58teHPnm1rj9ifvN35SkT3Qcw6w90sgv/JbyyvUZyEnJwcA4OPDOtVcvHgRUqkUXbt2VY6JjIxEjRo1cPr06YoUjSAIomqSdEx1XKy2RTDmP5On4gtX7PtvX+Xx1FNTdV7bLKgZnOycAAAvi19i4E7Vk/eT/Ce6LtNCYYXg8zMwV+6FCJ8I3nbNRFMBLgH6Jyp4wf77MMEMUlmXCotRkcvlmDhxItq0aYOGDdl41LS0NDg4OMDLy4szNjAwEGlpabzziMViiMUqjTI31/D+GEEQBAEgRW3bwK+O2aZVDy/UhZPICUEuQcrzl8VsOmjNapPq/gt8KDM4yrUtCxOPTDRGXA6afgjft/pex0gufWr1wdiYsfoHOZluvbFVKsyyMG7cOFy/fh0bNmwo1zyzZs2Cp6en8hUaGmomCQmCIKo4+yaX+dJ8Sb7Ovsd5jw1e72znjEb+jThtYw+NRdctKsvy4i6LDeZZcBA5AGCjHjS5lXmL95qUvBSsvrGat9bDtvvblMefNP0Eb0W8xemv7ladd86Z7WYi1MPA+lOni/7+SkSFKAvjx4/Hrl27cOTIEYSEhCjbg4KCIJFIkJ2dzRmfnp6OoKAg8DF58mTk5OQoXykpKZYUnSAIovIh01HuWVx2S2yxTHd9g9f/fd3g9U52Tlr5FU6knkBmcabynK+IlCbqzpF5kjyd4wJdApWRDAN3DsS8C/Ow8NJCrXHn084rj2t61FQez+swDwDg46Sdi8LFzsWgnABUmRyrgIXBosoCwzAYP348tm3bhv/++w81a9bk9MfGxsLe3h6HDx9Wtt25cwePHz9GXBx/GlJHR0d4eHhwXgRBEIQaFtgjP/1Uvx+ZvoUbYC0LhqwG+kIYFQS5qh4k+QpSKUgvTMcHBz8AABRI2Qi8c8/OccZsu7cNe5P2cmRU4G7vzsrE4xuhCOs0iGJrRqz/s6kMWNRnYdy4cVi3bh22b98Od3d3pR+Cp6cnnJ2d4enpiVGjRmHSpEnw8fGBh4cHJkyYgLi4OIqEIAiCKCsHvtNuSz5p8jRP85+iRF6CGh41DNZrOJB8QG+/MU/jxuQxUA+zNGSJ0AztFKiFikplUq0sk+ppnxXbHYqoi1NPTyn7qrlWMygnAFX1SUYOFGZaJGNmRWFRy8LSpUuRk5ODjh07olq1asrXxo2qdJ4LFizAa6+9hgEDBqB9+/YICgrC1q1bLSkWQRBE1cavrupYsbhm3DRpCjkjR/w/8ei9rTcKpAWcp24+BAZyNiisCpE+kbz9/s7+6BTaySjZFOGTJiVJAje0skimXZ+ijpfK6VOhLCgsCworBQBlVIdB1EtTz29giqg2h0UtC8bEvDo5OWHx4sVYvFi7ChlBEARRBhzcVccyMVubYM/nqrbI1wxOUay20H174ls8zX+qd/zyq8s552Oix8BR5IjfE38HoFImpsVNwzu73+GM9XL0wqGBh5RlqA1hJ7SDVC7Fprub8HGTj5Vzezp68jo+KpBDlY1Rc3shzCMM3k6qUtKKolGaZaoBwNFOO4kUL97hquOSIqAoC3DyKlMyLGtDtSEIgiCqEgenAlfWcdvO/cE9f93wnvv/Xfs/5fHhx4d1Rhoo0MyVMKTBEAyoNwDOds6IDYxVtvNZKLLF2UYrCoDKSvF/1/4PZ9POAmAfTvUpCpqoK0MA16oAAA5C1rLAZ71wFum3sijRrGY5OxyY7gXc3m2smDYDKQsEQRBViZMLtds0fRiM2Dtffm25wTFX37uKiU0nAgA6hHRQtu99Yy98nHzg5+yHQwMPYXk31VyGtjOMwU6gMoo/y2cTMelydlSv7SCA6oleU1nQVFYUCglfiudyl9HeMFiVsKmSQMoCQRDEq0R3/nTMZUEgEMC9dMtD3VyvnpvAw8GD4zjIpyz0rNnTpPtmibM4MgD86Z8BrmVAXVnQNV6BQiGQM3KOwqF+z3JRbLwVxBaosAyOBEEQhIXJ0L9VAACoHqu3u/fW3spQQ2NwtWcrVyqUBXuhvd7F1Nmeqyx81/I7rWJRplBUwjoqGqUsqMn1tIDrg6GZSVLderD0ytIyy4dGbwNXN2q3VzK/BbIsEARBVAWkxcASjZBzPkdGe90hjMk5yXic91iZilkfs9vNBqAKiVQoC3wVJNVR+AIoeDvybbjokckQM8+yJayLSitqam4n6Nqe0KxjoZnjQV1ZWHZlWZnlQ52uOjoql7JAlgWCIIiqAE8qY4S3A27v4rb584cuFktlyqd0Q1wcclEZWqhY6BXOhYZM9GYx4fOw4NICANDaMlC3ONzPvq88DnQJ5CSS0kwqVW6/BAU602SXrUKmtSBlgSAIwobIK5bi0ctCRAV7mLawlvCY4dVD9wAgsCFgr50jYNqOG1h1KhlCh3S41jZ8K4WiAKi2IRQYyuSozp/djC9NrY+ikiIcfHSQt+/H0z/ytjcPas5RHq69uMbpV3eiVGdtr7WmCacrykNXSm4bhbYhCIIgrIBMzmDW3ls4eDNd2cYwDKKnHcBrv53A2aRMPVfzkJ/OPY8dAVTjFm7SVaNg1alk2Llfg3OYcYu3XK56KlYkSDKFba9vw5IuSxAXzJ/W31RS8nTXCEp4ksDbvv72er1z6lLUnERGJmRSTqSjGufFVabNY2VIWSAIgrACO688xR9HH2LMmgsAgFvPcvHZpivK/qN3n5s24Yu73PMeswCPYKD1BFWbDmVBIMqHc8jfENoZdmwsfDwC+26kKc/LoizU8a6DdiHtTL5OF98c/8bosWKZGJMSJmm1x4fHa7WFe4RrtZn8fnVZhypZrgVSFgiCIKxAajbXP6DnouPYejlVeR7qbYLTn7QI2DqG26aIOghVc3p0C+C93CPwLG+7IoeCgvz7X0JWEIHklyqlQnNv38Oh4ov73cm6wzl/s96bOsfmSfJ4tyz4lIUAF+3Py3RfBh3KAt+2kQ1DygJBEIQVEKo9caqb9RUUSkyoe1CUzT2PG686Vndo9NHhkOB8T6vpyntXMCp6FFoHt1Y1MuyS4SBSLR2ai+eMtjOMEtmSTGk1RWffrLOzOOe9a/XG751/R9ca2lELfIqBycpCWGv+9vw0YO2bQPoN0+azEqQsEARBWAG19RbpecVa/Vee5CD8692Yt/+OVp8WmpEQ3X5QHTupPek/OgU+nKBdUVIRgvhWvbdUjQy7/75fbRtCc/FUz+RoDdzt3fU6hh54xK2OGewajA6hHXivMYuy4FMTaPURf9/9g8BSHcqEjUHKAkEQhBUQS1UhfueTs7T6d15hkwb9fuS+Vp8WmUncc6GaU52jmrIgzuW9XC7X4YQH4GHOQ+UxU7pkqMvr6+TLGW+p0Eh1FEWe+NjRfwcA4JuWxvsx6IIvIqIsPhpwrrylqRWQskAQBGEF8sWqbYaP11/WO7ZYKtM/WeYD3X3qoZLx3C0ChmHwxdEvUOzE9VmI8IpWHnP8ARh28awX6KZsEglFvHv7luSPrn+gkV8j3j43e1a2FkEtjJpLMy+DOmaxLACAnnsAYH1ObBxSFgiCIKzAsxztrQddPM8z4Az3Uo+yAABTMoHP7wPBTTjNqfmp2Je8T2v4t03YqpSFkhI8faGWnlnOPlWHaDhfDoocBACI8I7QL4eZiPKLwt+9/8bUuKlafU52rHIk0hWyqIE+ZYFvjjIpCxEGal/kZ5g+ZwVDygJBEEQFwjCsM2OB2HgHxqxCie7O3Z9pl6DWRCgC3Py1mmUMv8Xi9MMsfPT3RTSYsh+nLsaAkdtDktUCiiVDU57hUcOxuMtirIhfoV8OMzOg7gCdfSKhccpCk4AmOvvUC1YpKNM2RLVGwPiLuvu3jzN9zgqGMjgSBEFUELee5eKdP89AKADEJQZM02oUSvRsQ5z/P+557Aij531ZxF8DYs4+ta0HuTPy73CzIF5+nM05txPaoX1Ie6Pvay4EAgE6h3bGfyn/afXpysCoScfQjjr7FHUv1DHWYqGFXx3dfcnHgePzgRZjAEf3ss1vYciyQBAEUUF8vfUacoqkyCqU8ioA+ya2Q3xUoFa70WGUry3U8kvQRb4kH8P2DeO0ycT+KHw0yrh7lYMLyZmYvPUqcgqlkMsZXH6cBYkJypM66lUsN/fZrDw2ZFmwE9rh7Yi39Tpkvhf1nlabxRw4D08H9pffKdNSkGWBIAiigriSkq23PzLIA2M71sH+G9zUzQViAw6OADD2FBAYZXBYoaQEzvYizLswT6tP8jwessK6Oq/9pEtdLDqsnZPBVN5cdhoAIJcDmYUSHLyZji/iIzCuk56nbx2obwuo16ww5FtwcchFrQqVmtT3qW+yPHoJigbSrunuT1wHvP6bee9pJsiyQBAEYQPUDWC9+BuHemn16bQs5CnyHQgArxoG73HzaS4aTNmP6Ttv4vRT7ayNg5qH670+MkhlIn/4XFc1ReO5m5GnrI1hci2MUu5kqrZM7AUqxcHQdoEhRQFgK2rueWNPmeTiZUwCMPkJoKuMt7wEKHhhvvuZEVIWCIIgbIA2dVSJkVrW5Mbl67QsKJQFt0Cj9rrnH2QX1lWnkpEr1i5pHRvuptWmzosClWNj51+OIq+4fJUT1X0ffF0ddA/Uw63MW8pjoVC1pGkqAyFuIcrjet71jJ4/1D1UeWysH4RORHbs9yTX87lplhS3EUhZIAiCsAEUSZgAIMiTW9mwSFeehczShEnOXkbdI69YZaGQMdrWCgmjv5CUWEOO6GkH8N/tdB2jTUMkLL8vQHW36spjzdLZL4pUT+xjY8aaNK+gtL5D/7r9yyGdGvryLuz8xDz3MDOkLBAEQVQAipBJXdT2Vz3VT+5ZH83CvOHnxpqrecMsGQbYUhr54OKr3c/DnfQ85XGRTDubY4+aPTChs26/gSY1vLXaPlmfaNS9DcFXH6M8aFoWhkWpnDk7hXYyaa7t/bZjfOPx+KzZZ2aRzSC5zyrmPiZAygJBEEQFYChU8qNOqiJPQZ5O2DK2Nd5pzprAtSInXtwDnlxQnddoBWPILtS/beDh4IGPOupWFhpW164omWdCvgh9qFfcNIX9A/YDAD5qrF1/4Y26byiPwz3D8UPrHzCn/RyjczAoqOlZEx/EfKBlrSgzw3axRb2G/ssf6rplpHnuY0YoGoIgCKIC0JtYCYCLg/bPsasj28axLCSfBFb1Up3X6QZ00V1l0RikOU3QxH0gAMDZQYTa/q548Fx7S8LRrow5BozkwfN8joXFGILdgnFtGH+EwfTW07H13lYAgKPIEa/Veq3cMpqFmu2Ajy+xxzsmaPc/5i/4ZU0salk4duwY+vTpg+DgYAgEAvz777+cfoZhMGXKFFSrVg3Ozs7o2rUr7t0rf1gOQRCErbE0QZWSuV1d1pmxe4NAvBkbgohAdzQK0S6O5OrILs5H7mSoIiKurOcOih5YBmlUJn95iTuKnw6Ep12w2n1Ne448cjsDM/fcwvJjDyEuMSLMUwfLjz1Ekb4EVOXAQVg2B0qLk5PC325jUREWtSwUFBQgJiYGI0eOxBtvvKHVP2fOHPz6669YvXo1atasie+//x7x8fG4efMmnJyceGYkCIKonJxTCw1cOiQW/93OQOfIALjpWZgV9SNe5EvQYW4Czn/bFVBLQgQAqBdfBmlUC3LBg0kAhPi6Z6SyrXNkAK4+yYGzvQgyOQOJTP8WyohV5znnY9rX0jm2RM9cG86nQCQUYEb/aJ1jykpFF7symu4zgAPfarfPLd2WmpoNVEAlT0NYVFno2bMnevbkL6DBMAwWLlyI7777Dn379gUArFmzBoGBgfj333/xzjvvWFI0giCICsXXTfVk6+Zoh9djgvWMZqnp64pAZGKG/QqsLogHw3SBwE4jRt/ISAh1B0Khs5p/AGOHQ5PaI9xPtR//Ucc68Hd3RPu6/jh27zm+3XYdvaOrGXWfGXtuoXejagj2cubtl8r0OzJuvZRqVmVhTvs5SCtIQ31fMydYMhetxwN1uwEufsBcHiXrvx/Lvc1kDqzm4JiUlIS0tDR07dpV2ebp6YmWLVvi9OnT1hKLIAjCIjQPZ3MnONgZ/7M7sFkIZtivQFfRZfzP4WecT86CTN2c/sFxo+f6eIOqDLZr+FJVB2OHOgHcHA0OdkK82zIMoT4uGNyiBnZNaIsFbzc2+l7LjuqugmnISiE3EDViKj1r9sSIhsbXy7AK/hGAqy/QeIh23/FfKl4eHqymLKSlsclEAgO5edADAwOVfXyIxWLk5uZyXgRBELaOovbBuy0NZ1pUIBAIECp4rjwf8scxXH6oltcgsKHRc+26qiscT7+JWyAQoGF1T6WSM3uA4af+NacfYer267x9hmpAmFlXqFz0WwxE9NZutwH/hUoXOjlr1ix4enoqX6GhoYYvIgiCsDLFUnaRNDqioDgXWNQYEcInyqYGgke4/bh00e84GRDq/wl/llOEjNxivWP4HCv18XbzGpjR37CSsvr0I952qQHLAoNXWVsA8M7fwOcajv42UMLaaspCUFAQACA9nZv9Kz09XdnHx+TJk5GTk6N8paTo8CQlCIKwIfLFbI4DN0cjlYXjvwBZSZwmf0E2XAWli7+D/pj/YqkMcbP+Q4uZh/U6FS4bEmucPGoIDFgj9GHIsiCVMdieWLacC1UCgQBwC+DW+ri7j/035TyQ+5T/OgtjNWWhZs2aCAoKwuHDh5Vtubm5OHv2LOLi4nRe5+joCA8PD86LIAjC1tl0gbUQ8OVT4KVQ2/S83GE+egvPAABkdi56L3+eJ1Ye66taqcsR0VIYsiwAwCcbEi0viK0TP4t7/jQRWNEVmG8dR02LKgv5+flITExEYmIiANapMTExEY8fP4ZAIMDEiRPx008/YceOHbh27Rree+89BAcHo1+/fpYUiyAIwmo8zS4ybqCOzXsHAbvwi5N0J+5hGAbt5hxRnudLSuBTWqjJ01lVmdFP0Nw4WTR4vXEwavmXLZuhIQdHggdnbyD5hFVFsGjo5IULF9CpkyoH96RJkwAAw4YNw6pVq/Dll1+ioKAA77//PrKzs9G2bVvs27ePciwQBFFleaeFkX5WBjz9cuQu0LQtPMkqRKCHE5JecLMv5heXQFYaOrnu/aZ45wDb3tD5LeNk0cDN0Q7/fdYR9zPy0XX+UZOuNbQNQbCk5Yqh2JBnJIUQ+JbmXahu+raRObCostCxY0e9xVMEAgF++OEH/PDDD5YUgyAIwqoUq1VrDPAw8mGI0Z/JsMOVLrg7SHV+8VEWBixlrQ2hPs4Qud6Gg98RFD8ZilvPcpWVKw892aa85sP2DYx8B/yEeJu+hXE9Nadc93xVeJ6nUhYEMrEqIsKAr4qlqHTREARBELZGwp0MvU556nkHXA35LDz4D8hMAgpf6hySwXhBAntO2+//qTzoUzKL4FJjFexcHsGt3k+YuDERkhI5hALgfq6qjkKot5d+WQzgZC/C1D4N8L6ejI3qPHpZgO+33yjXPV8VEu5p+Kxkl0aXOJhWO8NcUCEpgiCIciCTMxi+kk133DzcR8thMF9cgoWHVAu5SKgnkuDxWeB//dnjUN2VJPMZJ9iLuPMcufNcx2gVPq4OSHii8mVwtiu/c+OINjUBAENbhXH8JDTJF5egw9wEo+b0c7PROg4VyHXHWDyQV0NtYWmobDqrZEmEzrDGp0OWBYIgiHKQkafKY/AiX6zVP3btReMnSzmrdnxGdexdkzOsEE4G0ybzoZk90l5or2Ok6YT6uGDWG9yETerb0Lef6U6g56ghl4ezaXI9ySrEPxef6A0RrWx0iApFF8k8HJU1Yhvu7AEAbLuRbRV5SFkgCIIoB6lZquiGzzZd0eo/rmlO1gejY7ETco3AdQTsloc+nzA+CqTaZafNyaAW3OyUyS8Llcf6aiH5uTni7DddsGoEG51RYqIi1Hb2EXy2+QrWnuFPBFUZYSuOChAkyOS0y+XWUYhIWSAIgigHqWqhkPcy8vEkS7VAZhdKOGMNm9d1LJL+EZxTJwGb4ElcGllw5HaGUfPkilVP922rtzUgS9mIq+WrPB64zLg6P/YiAQI9nJThnaZYCK4+yVYen3yg28+jsqGIGlkn68Jpf1N0zBrikLJAEARRHv4++5hzfj5Z9ST49T/XOH3f9TYQfaDLUtB7PqcOxAM5WwEyt5hVGpYmaBRuEpTwTlM3XFUfYkmXJfplKSOL3mmsPOZuy+g2LQyNCweg8ueQmWAx+b/jqiyXIhso5WwuFPkoVsvi8dKtnrLdXqA/SsZSkLJAEIT1kEmBc8uBnCeGx2pSlA3IrfPDqYBhGI5yAHAzJx68pUpnP7JNTSPKUutYJN0DgaH/Kk/HSicCAO6m5QMAzmnIYOd2m3eaZ/ZrlMcCCy2spvobAMB7cWEAALvSWhcyufHKgr1ItYyJRFVHWdh7TVVQcW3jdcrjlT6fWkMcUhYIgrAix+YBez4HFkRxn6rTrgEnf2WVCXXkcuDITODPTsDc2sCm9ypWXg1Ss4u0jAHqSYfUF72J3epCqC8SAuD3WRiwgv3XzR94Zz3w3nbcZdjETkNWnNUa3raOH0JqH9Zqryg0nRUV6HrrH3WsrVzwFZaFEpOUBdXEVcmycOK+ytdFXCLD1tbb8aV0DC768FSlrABIWSAIwjowDHDuT9V58nH234zbwLK2wMHvgWNzgR0TgCVxwMME4Adv4Ohs4OklQF4C3N4FSPVXVTQHvx6+h2Y/HURaDnuv9NxivMgX49azPABAdS9nDG3FPh1LSp3zNJ0PXeyNKCDFZ34PV/MtiOwF1OrIiWoolHC3HNaObommQdyoBAeREJ0jfVERaFosdl99hkcvC3AhOYt3vHooqZ1iG8IEB8cUNR8RzWiPyoy674dMzmDSfwXYJOuEHLF1HBwpzwJBENYh5SxQpGY+X92H/bdOV1Xb0dmq4zV9dcxzBqjV0eziKWAYBvMP3gUAfLvtGj7tVg+v/cbN0/88Twy70ifcLRdSEOLtjK/+ucoZYycyYiHTtKQAgFugVtOdH3ug5mQ2lG7zBdUWTsLnHQEAPk4+nPGrxvojvyQL503LzFxmOkb4I6E078O4dZf0jm2ltiiWxbJw8r7KqdHJXghxicz4MuA6EJfI8MH/LqJVLV982KF2ueYqKz5qzrDqn4dJ0TVmpOqoYQRBVC4elRZCstNIf3z/kOFr31kHRJSaY9f0Bab7AM/vmle+UtQdGNNyi7UUBYCNcjiXxCo+T3OK8eWWq4ZKO/CjmbUxpAVvzKFAIECvaDYZ8NQdqoyI4X6uYBgGtzJvccYvvDwHOZJs5XmUb1QZhDOezpEBBsdsfL8V1oxsgTZ1/JRtCoWrSFo2X5RTD14i8vt9+F85QihLZHL8vPc2Eu48x897+X0/KoIiieozWHEiSc/IioGUBaLSk5SThIlHJuJB9gPDgwnboEQCHJ7OHnf4Sv/Yzt8BHiHsccwgYMIlILI30GMm4FOaZpiRsb4PBkjLKVZGEBjL/hsqR7MbT/kTC33arR48nHQ79v3Ur6HOPg6BGtESw3fpHKq+yKoz9tBYXH3OtWrcyryFH06ravAs6LjAOHnKiIMRVpQGwR5oX8+f06a+JfHweb7J9334vAAMA3z/73WTr1XwyYZErDyZXObrzYXm9pKCDhqfWUVB2xCEzfKi6AWcRE5wM5ALffSB0cgozMDhx4ex9fWtqOtdt4IkJMrM1jGq4/B2QORrrP+BAo8Q4P0jgFvpE2r7L7Tn8A4HPjgGHP4ROPcHkHQUuHsAeHwaaPUR6xBYyujV53HolioXQcuaPlg2JBberoYT5/oYGPNwZi8IhQK0qOnDm844JtQLQ0r9GQwi0DCf2znqHPp6TDC+3aa9KJ58etLgbaq5VTNOnjJiyHcgJtQLbo7ay4+6g+KD5wWo5V/xdRB2X3tmeFAFoG5ZUOfTbvV42y0NWRYImyQxIxE9/+mJ9/a9pzdLnZyRI6NQtQi8seMNFJdY3uGNKAdyGXDzX9V5aHPgjT+BgauArx8DH50BJl5VKQr6cHQHes4Gwtqw5+sGAifmA/PqAHls2GJWgYSjKADA2aRMNPnxoNIXQRdHbmdge+JTAECv6CDUDeAuXqPb1lRGOIT58lcDDPEyof6CnP9pkg93DUvGjenxxt/HwhhSFraNbc0buqluWTAmruH/jj80VTSTkZvgP2FOrjzhr87ZMNijgiVhIWWBMIn/Hv+Hnv/0xLXn1wwPLiPP8p/ho8MfoVhWjHtZ9zhPSsk5yVh7cy3EMjaW/W6W9o/9okuLLCYbUU4ybqF4jlo2wmGl1gQHVyCqP+DkCQTUB4QmOKgJBEDfxdrte78EwzBo8uNBnZf+evgexCW698en71T5A/RoWA0HJ3VA8s+9cWVqdywe3BSfx3MzKzYO9dKa45aemghayNQyPg7fY/x1AFx5ntT5+KTpJybNWxZcHPR/f7pCSNVzJhiaAwB+2n3L4JjyUuubPZzcGdZEKDDSUdYS97bKXYlKyydHPsGT/CcYvGewyXnpjWXBpQXIk+Qpz8ceGosrz6/g1NNTmHN+Dmafn4055+agRF6C+9n3AQCh7qHK8WtvrbWIXNZCIpNotVnqs7c4VzfBqVitOmLNduaZ16emdlvhSy2LAh8R3+0DwzBIzy1GTqEUKZmFyC2WYtqOG5zaBkEeKkdMT2d79G5UDU4a4ZDf9KqvNX+vaCNM/tJiQFrEvgCgyVAgvI3By2LDvDnn97Puc84HRQ7Suqaxf2PD8pQTb5ey1UVUV3gcjQk1NTN8hcAAYOqOsvtAmJPyRnmUB/JZIIxCzsjR458enLaTT0+aNb88wzD4+9bf2Ju0FwAwKXYS5l+cDwAYsmcIZ+ymu5uw6e4m5XmTgCZIyUtRnpfIS2AnrPx/3t23dMezgmcY1XAUJsZOBAB8dOgjPCt4hrW91sLVnt/0batIHp9XltfNZZzxODUHDat7mmdy/0jguZr3evJxJNUyzknufkY+ui1Q5dz3crFHdqHKEVIoAGr6Gf6sm4d7a7VN7GrAh0ZaDCxpCZSIgeiBbJtmhIgOVo1ojl8P30OfmGCIZWL039Gf0+/l6KV1jYu9i1FzlwfPMmRxVFDL3xUPnxcYzOKoleJaA5mc0V8OnIdiHVEYT9SKhVUEurY+yholYg7IskBowTAMtt3bhnYb2uHj/z5GSl4KUvNS8ayA6/gz9tDYMj/hyuQyTD01Fd+e+BY54hwsTVyK1/99HbPPs3H19X3qY0TDEUbPl1aQhq2vb1WeP857rGd05UHxma+4vgIZhRmQyWU4nnoc97PvY8eDHVaWzkQkBRClnFKezioZjNn7bpe5il6BtABy9YyHr5V6+LurUip7PFEpAJs/jONcr6iQ+LXdOuSsGYIQwXMEgQ1dVCgKkYLHODFAhoNDg+DvYvipTiAQIPnn3kj+uTfmvxWDDe+3Mmw2fn4byEoG8p4Bp35l2+yNUxbcnezxbe8GaBTihSKp9oLm4aC9v13Ls5ZRc5eHap66/TR+7Ks/bFORmElfMSmGYTB7n/6wRvVMmsaiq9pluA5/FEthSp6JioKUBUKLaaenYcqpKcgWZ+NIyhH02toLR1KO8I795cIv2J+83+CcDMNgz8M9uP6CNefdeHkDW+9txY4HO9B2Q1ssubIEybnJyvFjY8YCAHb15w8diw2M5Zz3r9sfdb3rwknE/sjOODPDoEyVjS6bu2DKqSnKc3XHTlvm18P38NX60yhc1hUiRoYMxgs/xJ7GFnTFyZQraL2+LdbdWmd4IjXuZ91Hl81d8PXxr1WNYa2R/UUG5lRXhQUGPNkHAJjWpwGah/vgx75RcHO0w4XvumLWG9EYJDqMD+12oVlBAk44foIzThNQS/AUHshHstNg7HP8GiG7h6L2ps7Aj77AwalGy/hG0xBOwiGdaOZWAIy2LKhTwmg7Rwa4cJ1EwzzC4FSGuU3FWY+/wcBmoTr7AMDLmbU9vSzQ3n5TcPTuc519CvJMDJEFAKkOBWXHlaf47fA93r5rT3IwaWMinuWYz/qQXaT7vVsLUhYIDvez7mPrva1a7XMvzAUA1PWui71v7FW2r765Gp8f/RwPc7heySm5Kdh2bxtkpYV+jqcex1fHv8J7e99DWkEaLqZf5L1/lxpd8EfXP9CpRicA7I/bsq7L8E3Lb3BhyAW8HfE2psRNwZ/d/sTmPptx4p0T+CvmM/Te+Q1waJrSxJorMcGpzEbhs9qoWxP+79r/IUfM7zFtMSQFJg1PflGA+QfvIvrGXLhk3gQA3JcHY3znOmhd2w9OQdtQUJKHWedmAWB/eE89eIGLjzKRpWOxkMqk+ObENyiQFii3rADgdlouGv94CEsSZTgoY5VJ3/y7GCvageEHGwMzgjE0xgPXp8fDz40NSZxlv0Jr/ndFh/G+3W7+N3RyIVCYyd9XVvjmK4uywBNJ0Tq4Nefc2c6EyIxy0rKmD2+7rtoRCpxKFQ19loGDN9N19ilQZJA0BYkea8YvPJEzUpkcfX4/ga2XU/HZpism308Xg/48ozz+pItqG8sYp09LQcoCoaREXoIlV9iytW2C22DfgH0QacR9f9zkY4S4h2Bu+7mc9r7/9sXBR6zX+fm08+i1rRemnJqCg4/ZtlU3VgEApHIpum3ppvRFULCi+wpcfe8qFnZaiNbVuT9wbaq3waDIQXAUOeK7Vt9hYL2BcBA5INInEp4OHmj+7ycQZD0CTizAz4/uK99LZScz33C897WU4xUgSSkPE4CZwcA0T6MXzLNJLyGAHEPsVIWNdshbw8fVAQNiQwCB6unvQnIm+vx+AoOXn8WApafRfu4RPMspwtGUo9h4eyMAIKs4C7Frm3EyFE7bcQMZecXosVD1WXwvHQ4AiBE+xFf2G9hGaQEwp6ZK9hJ+ZaSa4CXG223X/aae3zHqvRtNIU/63nzTrUYvirjzzGw7E24Obvi6hcr64mJneX8FBQvebszbbqjapb0y5TN34X6WU6RUIIypAVHd23TFSNc2hC7+OKrym7j0mL/2RVlQ948Z0DREedyuLn8iroqg8nuAEWZBzsjx7YlvcfDRQdgJ7TCi4QhUd6uOhLcS0G4j67Ee4R2BNtVZD+348Hj4Ovvi2otrWHCRNftOSpiEVT1WYfPdzcp5b7y4AVc7V5xPO6/z3kPqD0GLai1UDSVi4OIq4Pz/AS/u4mWLL+DZ7SvYCRhAZA9kPwa8akAiY7B+8wYMU5vLtTRy4H72feSIc+DpWH7nOalMzgnpMpXHLwux4NBd1K/mjtFta0EoFKBEJsePu27ibno+pDI5UrIKIZMzeL99LTQWJaH5k9V46MQ+/fqWyPBShxf0ox0foO2Qw4BfXSD5BJCXxjrJCc38HJCZxK3NsPcrYMBy/rFZj4DkEzjn1AZf/XMD+x1Ui9U/srZIqck68bWu7QucVf04v7nsNGeavOIS9Fh4HPLwzwAADX1jMCthOxhwF5FVp+5j1alkTlsafCFjBBAJeH78FzUGJj8G/vtBuw9AL9E5/vel4NEJICAScNZ2ZiwTYh4nzLynJk9zO5O7hx8fzuZdUHdydLavOMtCWYs6KVI+S9UW7uupOXjttxNoFOKJHePb8iZ00iS3yHzbEApOP3iJuNqqraV5B1TWhmKp+Qo8hXirlLpQn4r7zvRBykIlI7s4G1vubcHLopf4MOZDeDh4mKUu/eLExdiTtAd2QjssiJmIlieWAje6w8uvHq61+QrymEEQ5jwBHp8FCjMhqNkOzYOaIzYwFidTT+JcGvsDO3zfcE4Uwqobq5RWhddqvYZdD1U+CCta/YgWvlFsyt5nV4DiXLaa4JFZQIlq/8/33FzgHNeSAXsXOEgLlYrCJXkdNBXeh6faf/blV5fj8+aGUwDr/VyO3Mfc/eyT5OXvu8HD2R4ioQDFUhm+/ucqGACzBzSCnVAABtw48UJJCYauOIdqKXsQKniOWZdfw6+H72N85zq4kJyFQ7e0Takz99xGstNgAMADdzfAzwdRnrUxoNgLSc+PYK2HB17YidBW6IET8lwk2duznvQCEZvyGAC2vQ98l6E3+x8vTy4C/9eZPR53Dnh4FIh5m8198Gtj7thrm9gFM/USILRjKyMK7dgsije2AQAC5IHY6+CACKGq0NHdVnOxpDNrVvVzc0SAqztelrBP0QKH52Ak/hAIVMUXc4rz4V567Zhty5DnqF16WSAqAiNzV55f+r4b5AyDSTPHYpHDEu33Kc4BUs4Bp35TtY07B1xYCZxdqmqrGw/EfaRdwOq/n4BrW4Bx2uWhy4ScZ1FrPka7TQ+F0kJMPz2d0+YgYvf+hQK13AUVaFkoq4JtV6roqkdDvFNqlr9amqgoRM1q0CjEEzP7R8PDyR7t56p8qyZvu4aexoStqqFvGwIABi0/g+Sf+UtEmxh4oRd/d/b/rpujHef3ff8Nw9svloKUBQNkFkhw6FY6fF0d0KaOn1ZcdUVyO/M2Ru4biTwpm4Pg8OPDyBZnI9g1GN+1+g7NgpqZPOfV51dx6PEhrLy+EgAwNep9dNw5mTXZAsCLu8D2cRBuH6d9cUQvCBu9hZEOb+CB4xO8FLNPQ7q2AJo/fIIclzo4Xngf39Ybg+ab3gdk/HHNJQ4eEEa/CeHFv/gFlxZyTq/Ja2KU5HMscfoNzvIcFAmFOPVgN1AOZeFZTpFSUQCgM7lPw2BP/O/MI9gJBdg+vg3+PvsYP++9jepezpBkP8M/Tuyi9ITxw05xa/y8lzWhCwUCtKnjh3xxCTrWC8DdjDxkPUwESj++h/Zs+Nmlx97YnfE2gHeAdDkcHMSI6lOIEzfnIl0k4ioKCn4OA8Lb4nlAG5S0/FCvd7oShaIAAItLLT17edIsKzis9mSunpGxlHAh94ct8919mFyXW/vA28kbL0sfrIV2uZBJ/LFrQlv4uTmi5czDENipnrrVFQUPu0CUMIUolOVhxpu1MOvfLNQP9sCfQ2PhVRrjv/CnGZBvToUw+Rgw9F/W6rL+bXaCFd24wnrXBHr+zFUWHFzZapajDrLKxYFvVX3P9XvimwRfpUkn07L06fNdCXZTRYe42Vdc+mR13wShADDWwV9hWdh66QmGtQ5H0osC5ItVvyn3M/KQnsv+bsRHBeKPofy/e+qhr8YiNWIb4t/LqVh39jG+6slNyGXOtUFh4ejXJJjTbk6FxFRIWeDh8s1DcJOk4/GzF/jsvDdyJCLURwq+tQ9GTJ0a6Ox9B/5ZB/FEFokbzm/Czk6El/kS1Atyx8DYEIT6mEd7f5T7CPmSfKy7vQ4X0i4gvTAdMkYGZ6EDiuQSZVjdg5wHGLF/BJoHtsKoiK/haucLL2d7eLnYI6XgLnbdPYmM1MZ4kcea5kJ8XBAV7AGp81msuD0bDNj/IM3ya6DnzqkAI0Ya44105zqIlt2EUKrDqe3OHuDOHrQBsJiphqEhfpA6sD9aG7zaIbs4AydCGmPbvZNol/Uc/XN24DUAV5wcEZv0PW86V8a7Jm4W+WBCziBUz2iENJd4fJs3Ex1FV/BrST+0Ft5ErsgHZ0WxGCFdB1dhCVz9QvDm2wuw7n9J+CB9IjanTMKbYd5IKn4OacYt2AfUB8MwePC8AEGeTkaZMAHgO7W8+0LIcdDhC9QWsp/5VOkwrJZ1ByDAjD2q/fPoaQeUx6nZRZgg+k95/pvD7wiS52KYYA/c7OQQeYfB3bMW0P9PVmna8wVQ8jcuODliuacHMkq3HjKlqj1LQAiJxBn/XkoFnICCWh2QOXQxfGQv2AXmv5+Ac3+ylpn7B+F//yAen14OZtifEGQ+BIKbosinPn7afRM9GgahXd3S+gmZRqbNffMvYMvIUlHsgOCmrFUoKBpw9YcUdth36zn6iM5wr+v2A3zqckMXZXIZ0iUqZUwgKsShSe1RJ4C1EkzrG4EfD/BH2gS4uaNQKkJhQR6iQ5xxbXoLrTECoRCCt9dwGzVzMQDAJ1cBu9LsDx8cB1b2ZDNItp7AtoW2YF+txgL/689aTwDWr8CYlNSG4HMarda4XFN+3kylJKuHTxqqs2JO1C0LpkQCCkufpK88ycGLfLFW9sSu81XhsPqsF13ra5f2NoS+cE0FEzcmAgCG/B93u8qQ46YpKJQFO43txLK8J3NhE8rC4sWLMXfuXKSlpSEmJga//fYbWrTQ/s9fETAMg4/OTkS+QoULA9wAKNL9nCoBTimdbO/ALXcn8qTVIJM74XRhDtam5MFZ5ABnkTPkkMFB4IS63p3QvWYnhPo6IDU/FZfSL6FXrV6I8mmMpBcFCPN1gYNIiOISOSQlcrwozMKfNxbgwKM9yoVcQT2JFMufPcETezs8sLdHhESCte4+2OXuiPPpZ3Du6SAUJI0DIAAj9YFLzUUQOaWBkTlB/LwXGJkTbhU/x3/ZSbBzZZ0BnQqqo2/hC3yTfwJCABfk9fCpdCxSxIEQQo5m7tlYPigSnt4BgFcNMAUvIPtnDOySVCa/KMEzHEhLx3j/ENQrliMq6W8AQJtbB6FeU9ABQPNi7n/+BFkMvpKOQTp8ULvEFQ+y2R/Ph6V124fjK2wY0QoLl5/BfDWjxS6vrjjwaXsIHe3gCmB4GwaTt+bhzcL5cJL/iGKhADdWvIY8/zewNKc1Tj13gKuDCK81CkZ2kQSFEhmWDonlVR4uPsrCvTvX4AFX9GpSEx/cGo6aUDkcTrdfjS8c/sEbxd/jLqMeCsZAkdV+iOggPrPfwpn3W2Hp4iUD8CITeHEZCGoEHJsHlGatHFGN+4Mgk7HOToveaYxWtXzRcW4CkjJkcKkBnEl+iqYJ/6F9PX/EhBSimu8ERI/8CrXyLiJn13cILrqLGkwqsEplOnUG0FnWBCcuRKLdO13ZFMP/jtX6DLRo/C7QcADQoD9QkAG4BXLKJz/PE6PlzEOQM8Bn0rG46zQMDASQfHQBjgF1tKa7l32Pk6lz1pu1lYrCr5d+xfK7y+Hm3RV8P98igUgZ+VJQYkKEhjOPh763WpGnao2Ab1L5rxWKgDdXAnNL8xTMq8umqy5vFspzf2i3mbi1KNXYyhgWpfLkqcgICHVEQgEc7IQm5ztQf+dZBRK9uVw0lYUG1TxwszS9toeT6cubKbLaaTzmm5oASh8KC4fC72P661FYdSoZU/o00HeZRbG6srBx40ZMmjQJy5YtQ8uWLbFw4ULEx8fjzp07CAgwg9ZuIs/S0+AiZ8CUulHJIIBMAMjU/vN6yuRwlNkjx64E+XZyCOxSYQdAXvqSQoxcmepHMPX5WiQ856Yg3nR3EwRZvZH7vCmEclcwUO3TeocuQ4lbMmd8+8IivJubh1ZFxUhjfNBInolGYtaZb+bLdLyT74Ch1QIhFxXBrc48AICsKBQiJ7a8rkBUDKdq3JBIAQO8m1OAL7NOQwCgROSME/W+xMe3oxBVwxPMy0I8ySrCuTwfbHrii6GhIcgvkKDZT+cAjAEwBvYowa+Oy9BTcAp+Mjk2pPEnQ2IgQOFbG3Dt3FG0fLQMxR2n4pBbX0RW94FdrgTpf7Fa+oPn7A//iDbhyjKx7ev5o1UtX5z/tiu++ueqMoXvj/2iOOlh324WipP3X2DX1WcQiQMB53S8EBWg65M/0Q5/osDREU3Ff2DjBVWmxw6zD+Pcd921/qP/+sdSHHNkE0RBR/p5NyYff9a/giUucZjQuS7un96BVucm4B+/DzHQ7gQc0y/xX+jix5Zd3v0Zu31wSH/sPiNzwraPWqNJDdahLj4qEDvvsF7YIudUQFiIY3ef40zmOohcklGUMgxgHABMxZd2G/GRHRtumQdnuIP1BekiuowuosvAP+u5NwtvB8TPZJ++XXyAkBZAURaQn84qNQDrPOkepCXnxI2XlU+QEtgj/7PHcBNK4ejK78GdVpDGOc+RsFapopIiLL/GOk/KPQ/xXlsiL4GrA6tEFWpsSenl8SnueQT//rNOXDXyJqzpC4w/D/jWNm0eM5NZrDs6RZF7BACOphzFl82/rAiRAACJU7qhRM6gkZrFzRDqe/Qbz6egWwPdT9OaC/b28W3w2+F7+PW/+xAbYSXQxJDPgjpyDSWmUEeVyLLwsPR3UBGdMax1OIa1Djfb/GXB6srC/PnzMWbMGIwYwWbrW7ZsGXbv3o2//voLX3/9tYGrzU9wUDXseDcRa/+7CFd5Pt5N/RECkT3Q6iPIQ5pD/iwRdiHNUezkD9zcjPP7JkIqAAoFAnjJASeXMORKpZDIiiCACHnCXBxyscctRwc4MAxeikSQlv5nYLx3o7rbQQjya4FxfgI7iRfyiupA6pYMAcNgRVoGfGWlf4BiPzxgorDILQ7rpR3gZS/FoNolaFi/AZpLzyP6wirMz7iKiYGqsrwiZ3ZRdJAz6FFQgOd2IhQJhAiVliBCIkHnwkKElsgAF18gKBp2veahvV9dJJZezzAM5h+8i9/+u4+dV59i9elkrbSnUtghu/cfQJQLkDALuLYFcv8IXGwyCzN334C37CXejfVHl2aN4FqtEVo16AFgFpwB9Cmdo24QcGVKd3y//TrOJ2eiTR0/TO0ThfpBHniWU4x3WrBP7r5ujvi6Z30kvyxEv8bB6BzJ/RERCgX4fXBTTOyah2E7tyEX6ThcvTfqPL2H8IIrcBWIccNtHDaFTUONpE1oKz8PyIGXM/zhExQOgb0TYOeEc8mZWO3AnwcCXyUDds7Ahb+A/ZMRnnEYcyYtBq5tRuj5DwABMOTlr9xr3IKAT28Am4cBt3ezRY8ierCVEle/xi7ECt5cCVzkKg//fNAFjQNUnvcz+kej2Y1izC6t5dW+1SnUd3gP69LZPX0791soyY0BIMCckncwp+RtCMFADiEaCJKxweEneAjYBfYiE4mmDikQCISAVyjQbwngVYN9wlbg5MF9+tbByfuq5EJ1A9zg5q4/EuVlETcZUXZxNgDgyGP+BGDq1PKqpawumpqvwxLAR0AUkFFaHKpBP+D13/QO5yWsDfCotLgZIwPW9AM+LUdhNc8aQE75Mo5qhk2q46jm6FoRqZ7VcXFglxh7kcAofwCAa1SxN2Da1+y3FwkR6MkqR9IyZHDUzM0wtFUYJveKxO//3ceShAfwcLJDbjFr3pRpKAtFUhkYhjGLw7nC+fmvk0lWtSaoY1VlQSKR4OLFi5g8ebKyTSgUomvXrjh9+jTvNWKxGGKxyoydm2v+5DuuTvb4oFer0rOuKtkACL3YhcsJAGLeQbuaHdi/7pxU9unC2Ys7WWEm3rzwF8TJZ1EMB0Cai9vPz+I3by9ccXJErr0E8C7dQ3XKBzxYz/GOxSXw7LIO9wrsIbR3RO16DdHc1QHdXBzwqZbEdSGIGYQuZ//A4qNT8djeHt4yGb7x94VcIMBv6c/RulgMqG9puAUBdToAjd4G6vfhNXsKBAIMaBqC3/67r/RCVkcoAOa/1Rj9mlRnG3r/AvSaByGA5gIBtjWN1bpGF54u9vh1UBNO21vNtTO91Qlww6FJHfTOVSfAHe/GNsXSK2fhWDMU4e+uANa9A9zdC7uSAgx+wHXa85U9B1JVPxI6N8BG7leFyzUfzVoE8tOBoz8DR2drjw+KZj3qm40ARHbAW2vY8R6lTkv+9YCJ14DHZ1g/AGdvMHW7aykL7g7unHNXRzt0qBOiVBYeFV3C0JheWFeqc3zRKwDDonrgeZ4YdiIBtl1OxZx9rG/ATSYcjcT/B1cUoQiOkEOI4U3DMK1vQ72fqS5e5ItxITkLtfy56XA1v0s+NJ+Gs8XZAIDnRYaT6XzT8hvsS9qH46nHMef8HAS5BqFbWDeD12HIP8D8SPa493yTHQkBAJ4af5c5j1knRVEZ6yEoFAWFItPH9Kqp+pQFdctC97DuJs9tDhztRJDKjMt9ov5LFBnkrnMcoMrJwGkr3ZowxUqgYNtlruLp5mQHFwc7DG5ZA0sSHigVBQDQzFLOMGz4pL7slZUZqyoLL168gEwmQ2Ag9wkxMDAQt2/zexvPmjUL06dP5+2zCh6loTk8ZlkArCm3/edwbA8o9PuWAFo+TUTW2SX4+8UFvJQVIcLRHy+keTiGQjyzE2FMozGo17wr6hkrh0AAtPoQ7cNas3nmvUJR7cAXeFaUidbvn2HlkxSw+9MuvkaH1YX7uWJoqzD878wjZdu7LWtg+utRyCmSwtdNYx4zaNXmINwjHABw9cVVVtsfvAHIfw5sGQEkHwcCowHfWtiTXw/r74vgBAkcIYUjpHASSOCDXAjqdcck8TI2nLPvEqBGK9UN7BzYkM/nt/kVBXtXYMg2wE1l6YFQpFIUlPM4ArU6AJ9cAQRC5DLaHtx8HuzeTipLQ2ZxJiYlTFKeP8p7CCd7kdLR9qOOddCmth/Wn3uM0w9fYnLP+ujRMAibLqTgyy1Xser0I7Sq7YseDU0LMwOAAUtP4dHLQs42zvoxrVC/muFFWKEseDp6Ikeco/ToN7St8GunX+Hn7IduYd2UtUQmJUzChSEX4Cgy8HftUQ34/iX7XZT1b9WD53MqfKn7N0Af6k+nUf2Ad/4GvMNNnkafsiAQCLAyfiWOpR7D8KjhJs9tDn7sF4VPN17BuE6Gt2seZ6q+f3GJHPrsEXwRCApHQ0M5E9RhGAYpmdrpmgc1r1E6p/Z9+JSRQkmJWZWFrvUrfiteF1bfhjCVyZMnY9Ik1Q9jbm4uQkP15xq3SYIbw7v/nxiv0Ty+OAcoeFH2PdBqjZQm5KbDNfZ7y5iQ5Ye+UegTE4xa/q7KNLkAtBUFG6JN9TawF9rjXtY97E/ejx41e7AL93vb2QXePxIQitBJIoPsVjruZeTjzIOX8PdwxL3MQtQLdMf016MA9AZSzgK1OmnfxK+eyrM+KJr1pM+4ySaTqttdqSjMPDsTD7Mf4vcuv+vOy+/IKgRpmdrZATUtCwDr1b65z2ZsuL0Bux/uRrGsWNmnKNutTkyoF2JCvThtA2NDMH3HDRRIZPhw7SWcmdwFQZ7GpxlefOQ+HpWWcFbExA9vHc5JWqOPl8XsNkRNj5pIfJ6IhCcJeP/A+wj3DNd7nUJRCnQNRJvgNjj5lN0SGH94PJZ315EoSh1ROX/2PKprt20cCozmD63VC6Ox4PCV2jYCfcoCADQLalam0Gpz0b9JCDrUC4CPq+HS1WeTVBanL7dcxcb3W+kcy6csOJRaFq6nGm91Hrz8LE4/VG2LNanhhRXDmivldbI3LtKhUCKDcX/9+nGyF6JYKsfErkY/LlocqyoLfn5+EIlESE/nxmOnp6cjKIhfS3d0dISjo+0uUuXGyZN92RACgQAtdOR5t1U8HT0xJnoMllxZglnnZqFltZbsIiMUAYGqqnfODiL0iSl92ue1YrsBdbrw36TtRAAMGxXQfDT7pBoYpap+CEAik2D9bdaJ8OCjg+hTuw//XKWcfqq9/abLmz3SJxLTWk/DuMbjMClhEpJyk5AjzuEU5NKHQCDA4JY1sPx4EgBg/sE7mPNmjLI/q0CCg7fS0TDYE+l5xegUoXrKySmScnJQAIC7ox0m94o06t4AlHUdwj3Dkfg8EQBw+tlpnE/Xne0TAGp41FAez+kwB23Ws1lFzzw7o+sS89JwALD/W26OkCcGsj7qQq7mFOdWtrA4hmGw7f425fmgyEFlk8XCGKMo8CHUE2XAZz24ncY6l+eYkMFRXVEAgC6RXMWGz7KgoJafK7IKJcgqlJrNydHJXoRiqdxoJaUisKokDg4OiI2NxeHDqmQrcrkchw8fRlxcnJ4rCcIwo6NHo45XHWQWZ2LO+Tnmv0H1WODttayvRkB93iEpearIi/RCVikukZfwJq4Sy8S4/vI6p+2v+L8MOkz5u/jjf73+h9392eJHRSVFWqF0uviqRyR+KC0ZfOoB9wdz/PpL+HLLVfT69ThGrDyPC8nsE196bjGO3+P6FdQNcMPid5vq/VFV55la3Ytqrlyzvr66Hh80+gA+TirF1cPBA1PjWB+PMA/DTphmwcUH+D6D9S1oNqp8c6kn02rQV/c4PTRa00hZqrtrja74NFbbq6kyI9OTpIGv68FznvTZJqKZqtpepPv/oJ+7ozIqq1BS/po0WQUSZUIpoY1s6wI2sA0xadIkDBs2DM2aNUOLFi2wcOFCFBQUKKMjCKKs2Ivs8UPrHzBk7xDsergLQ+oPQZRflOELzUhyTrLyOKMwA7mSXAzePRgMw6BfnX4okZfgw5gPUcKUYNDuQbiXxZbB/b7V93gr4i2T7qXu6f4s/xnnCVwXdiIhOtYLAHADL/LFSm/u0w9ecqIbADY/f2yYNzrNS+A8QelKf6uPJ/mqFND1ffgVLQWxgbEYUHcAEjMS8WHMh1r9db3Z9NFSvkyIliR2OCDOAy6UVq78pT7w8SXTtvtK1KwTZqjb8EHMB1bLq2AuNLM96lMW+HIwBHup3n9GbjECPEyv4Omgkb9Bn8LuZC9SVoM0h2VhkVopbHPmbigvVrdxvP3225g3bx6mTJmCxo0bIzExEfv27dNyeiSIshDtH40e4T0AAKtvrK7w+yflJimP9yfvx89nf8aj3Ed4nPcYv17+FUuuLMHjvMfYl7RPqSgAKFMBLHuhyhtfs6qnPvzcWXNrsVSO3OIS7LuehkHLtU3603bexJrTj5Q/iHUC3HDk844mywmwfhwKmgY21Tt2VY9V6FO7D76P+55Td0SBuz3r06FIg16hOKr5k+Q9BU79btx1j04Bx+cDEmWua9NrefAQ4R1heJCNc3kKN2JDv2VBu+9ttSiqo3cNR9YUiLWtAY4mpG52thfC2UFhWSi/svAiX6VAkrKgwfjx4/Ho0SOIxWKcPXsWLVu2tLZIRBWidy32ydcce9rn084jKSfJ8MBS1C0LmcWZ2Plwp9aYJ3lPcDfrLqeNz6nRFDRzGOhDEQsPADHTD+DDtTpyTACYuuOG8nj3x205pXSNhWEYpRNmA98GOt9rsGswZradydunjiKFcYG0QG+2P4vRW00xU6SCNsTKnsDh6cCV0qRYuhxfTSDSJ9IsMf7WxtOZG4KqT1ngc7JWd3rccvEJFh66q/Pv4sHzfAxYekqrXdOyoA9nexHulvpJjFlzodx/g+pbIKQsEEQF0iSAjfnPEmdxvMYzCjMw4b8JOPfMsHMawzCIXh2NkftH4vV/X8eD7AcGrwHY+h6aiAQiHBhwAC2C2IwO2x9sV1bmBIAg1yA09Ctb3gPF/n+ou2kRQs56nqR6N9IOFfyud32j/RM0KVKrKBrgEsCpiqjO7PazDTqEAqraB3JGjozCjDLJVC6ajwK6/cgem7qVkH6T/VdUNuc/db5v9X2557BFSvQoC8PiwrXa1HMvnE3KxMJD97QcGLMLJSiRydHll6NKh0h1TCmv7WgnQpFUZVHI57FUmIL6/yuRDSl/pCwQVR71J9dOm9gQyEJpIV7b9hoSUhIw6oBhJ7WbmTc55/229zP49J5ZnKn08lcnxD0E1dyqKffaFVEBAPBZ7Gc4MOAAp/iPKcxoOwMAkJCSoMxwaAwbP+APT1v0TmP8zpNgqW5g2S0f+VKVA9qXzfhTD0d4R6CRfyPePk2c7JwQ5cv6oijCKCuc0mRtvEWhNMlXM43L2JTt5dmCcLVnrTtejl5lnsPW+LGfSlm+k6Y7BJIvpwFfcakX+RLl8cGb6Wj8w0F8v/261jgFphSFeprDzc9gai0MTdSdKfVFglQ0pCwQrxxyRo73D77PecKVyfXvNW66s0mrbcbZGZzze1n38N7e95Thjx02qjJNft1Clbo8xI2tIlnHS7u4kquDa7lMybGBsQh2DUaeNA87Huww+rpGIV4IcOcuWP991gGvxwRDIBAgujrXh8LYyp2a3Hh5A102s6Go7g7uCPXQtoB80vQTrOm5RqfFgY8OIexnfTTlKG6+vKmMDqgwFNUcJUZ44t/ZrTq+vYv9N+8Z/1gDMAyjdOwUCatO5sDGIV7K4xP39eeQ0MSOJ3JBfWtgxm5W8V9/LkVrnAJTLAuauR70WUIMIS6RYc1plTXSFDksje1IQhAW5NdOqnoNSxKX4MrzK5z+C+kXOOdimRjvH3gff179EwA41REVHHx0kLOFMe30NFzOuIz3D76Px7ncXP/qse8SOfuU06d2H+Uip8DVznQfAHWEAiHeqPsGAODHMz9i9jme7JI6UM+R4OfmiFr+bkrFpX8TVSIikVCARiFlywWSkJKgPFb/TBVWlgF1B2B09GiTaxjE+LP5If5L+Q9v73obK66tKJN8Zaa0qBXE+ezWQprup1Zc2Wi2274sfgmJXAKhQAg/Z/6CXZUR9b36IhOdBvWVrQb0Rzbom+O73vxRO5o+CuWxLGy5+IRzborvhKWxHUkIwoJ0qtFJGRWhUADUUV/EAFYROP3sNH67/BtkchkOPmKz833U+CPU81ZlVbuUoaosqR7N8OUxrnldKBAqU1Ar8vM7ihzxa+dfMbmFqjaKoeyFxjCg3gDl8ZEUw0WZFPRvEoL1Y1rhjabVsXY0t0LGkFZhEAjYH/HTkzsb/EHWBZ/SBQB/dP0DXzb/EpOaTeLtN4SvMzdv3q+Xf9Ux0kIolJvMB8DSOGBZG6BYu54Kkk9qV74sB4rU2M52zoZTXVci1E3xV3jq0gC6F2/NSpSAqnojwK09YQpvxobAz80Rb8aGcNoP3crAe3GqHB9ltSysPJmEb7dxlUy+DJXWgpQF4pWhSxhr/mZ4ss2vvbUW+5P3K8/V9/sb/6+x8jjKNwr/vP4PRkSxeUDUFz/1bY0bL1VRAwpW9liJue3ncvInCAVCZbSGvdDeLKFvfs5+2N53OwAgV2JaobW42r6Y/1ZjRAZxfSYc7IQ48VVn7Pm4HQLcy+65nyvml8ffxR9DGwwts6+GPU8RJ75smBaDbxvrhXbabaTrsTiUAUWa76qkKABALX/teijqBHs6YXS7Wrx9fBEEnP/xRmgLfAENXi4OOPtNF8wbGKPVN+U1VWXIkjIUsAKA6Tu5flFDWhnOk1KRkLJAvDK0q94ODkLW69zV3hVbX9+KDb03KPs/P/q5MrOgrifgWp7sD5Qii+DT/KcAuIqCJnYCdn/fz9kPPWr20NqL93T0xN439mLfgH1m23f2cGQX3TxJHp7kPTEw2jiqezkjwkAVQEOoKy/jG2tWRik7ziLtKITppyuw4JwLTzr07GTjrw9vV6bbSkodJKuasmAoZPCLHrqVaoFAgPb1/DltxjgsRqg57fI9UKjL1bSGl7KtVS0f2ImESp8fY0txGyK/uPzZIM0JKQvEK4OrvStaB7cGALQMaom63nVR37c+ApxVNQ+eFbDKgq4QvBD3UudEb9Y58XEe65uw8OJC3vEtg1pidU/DyaBC3EMQ4GK+CnPejqqqlOfSyli3wAIo6lb80uEXfBDzgdnmreamHd6pnqTK4vAVgLp3SLtN12OtoiCZiSgsYFVNWTBE/yYhevs1czWo15DIKpBoDgcAtK+n8vkwZD37Qa2ke9Ma7P81xdZciWbt6jIiEtrW8mxb0hCEhXm/0fuI9ovGyOiRANhtgPWvrVf2KyohbrizQevat+qptg8Ui3GOOAcMw2Dd7XVa47uHdcf/xf+f0SGA5kQkFCkdHVPzU5EjzsHoA6Ox7d42A1daDqlcqsw7oXBINCctg7jJ3Co8IiLqDe75lXWAZgpqnu0SAECB4UyDfIhLi1nprGb6iqLpt6CuLMSGeWsOBwDI5MDKEc0xs3+0QQtaVLAHgkrTSI/vzD44OJYWfSpLFke+RE425NsIwAZqQxBERRLtH411vbkLe4BLALrW6IpDjw8hozADF9MvahUzcrFz4YQ/KtIx54hzcDuT/6nQHM6K5UFRWGn51eXIk+Th7LOzOPvsLPrX7W8VebKLswGwCpq/i7/+wWVAsfWioMKVtAZ9gRtbuW0v7gGBqv1smDm8UeGz4GCGpE5VCc1tDPUIBV9XfiuMnGE4lVX1IRAIcOYbbjVaxXp/PTUHrWqZVqh60qYrWm22lL0RIMsCQQCAMuzs53M/Y/i+4Vr9HUI6cJzoFMpCsaxYmVRJM29CTU8e03QFoojaYMDw5omoaLLEWQDY5EGm5FAwFk3nSEO5M8xOg75A/Cyu/8HltdwxMn4TOAZvLtMtL6dfBqDfZ6ayUidAv5OjPjQtC99vv6HcfpDpSMesL620MSS9YBNy/bT7lsnXbrucqtVma6m7SVkgCACBrvoLlyme4BS42btBJGCfEjfeYePm+9bmlhiu6WFdZaGxf2PlsUytFHJaQZoVpFFZFtT9KcyJwiSvQPM7szgCARD3ETB8l6rtzGLVMcMAuz/jv7Zed/52A6y+yfrDOImq3jZE+7pc69M/Y1l/oyZqzoW64EvMtOpUMgDg9AP+zKu6lAhj8SutUxHsafx3cSUlG4sO3TM80AYgZYEgAPSv0x+9avbScjKcFDsJIoEIo6NHc9oFAoHSulBYUgiRQIRetXpxxii2AayFm4Mb4sPjtdrX3dL2r6gIbmWyT1xeTl4WmV+RC0OBKemuKwRNS4dzaQRFGa0K6iW5+9XpV0ahbBf1XAv+7o6IDfPGzR/iseXD1gav5atbkl3IWhZSs1krjGZisRGtw8shLfB+e/bhwNgtiCKJDH0Xn8SCQ3d5+8ubNtrckM8CQYBN6jO7/WwwDINGa1R73cOjhuO9Bu/xhjR6OnoiszhTeRzgEoAA5wBkFLGRFIpqiNZEPVGUAvUyz3JGjjPPzqCuV12L+BEoOPX0FOZdmAfAcpYFR5Ejx7pQ4ZYFdWp3Bh78xx4zDGt10LB8oHZn4M2yZ5pML0xXHg+sN7DM89gq6taBYaVJj9QrpOqDL1Qyu0jKcSSc2T8a9QLdYS8SoEAiK3MKc6W8QkU0hHEWimP39Du1FksreBvNAGRZIAg1NPcJBQKBztwHng6qJxOF5/2ybsvg7eiNd+u/azkhTeDL5qpMkoptCUXEBwDsfrgbHxz8ABP+m2BRORZfVpnjL6brLoFdHsY3YfM2+DuzSs/ljMucp+8KpXqs6ni6F/DwKFCioSx4Vkd5uPr8KgA2Vbat7W+bAzu10EFTayTwZT7cnvgUX/1zFR5OrFLg7CCCg50QAoGg3IoCoFJujPV9MFRRUmhj3ykpCwShwfCo4UaNU0/coigNXde7LhLeTuBETliT5kHNUd2tOur71EffOqxPhfoTqSJr5Y2XN3jDt8yF+hN/bGCsnpFl552Id/Bv338xs91MZduWe1sMXnfsyTFsuK0dKguwIW0F0gLTnSU7fsM9X/sGkKvmxBYzGGinw3/BSK69uAZAO2S0qqCuIJhaI0FX+OKmC0+UizlfWujyoIhekBqZwdFQGgUbC4agbQiC0GR8k/EoLilGpxqd9I6L8I5QFqQa13icst0Snv5lxUHkgN39d0PGyHA+7TwAbsKpEkYVIporyVX6YZgb9VBUSylSAoEAtb1qcyIDUvJ0VxYEWCVm3GH2u4v2j1aWugaAAmkBPvnvE5xNOwuAtVgs6rQI0f7RhoXRXAnkJcCB71Xn/ZcankMHckaOA48OKMtx1/aqXea5bBn1xdzRxBoJFx9l6uyTKpQFMycysDdxG8KQ5cDWLAukLBCEBo4iR3zb6luD48Y2HosGvg0QHx5vE/4JuhAJRRBBpHTeVFcW1J0A0wvTLaYs+Lv440HOAwAwa6ZKPtSTMRnKbJiap3raP/L4CByFjqjjXQeF0kK0WteKM/Z50XMM3jMYV9+7apzZ/6MzwBK1OR4aX9RLH3uS9mDycbXiY6UFyqoa6oul2MT9+3ea18D55CzePoXjoDEpoE1BYVkwVlmQG7Lk2ZauQNsQBFFW/Jz9MKDeAJtWFNRRLNI54hwUlxSjQFrA8R9IL0jXdanJaG5pnHl2BgAwo+0Mi++vq5dqlujKa1CKQoEBgD+u/oH+O/ojR5yDY0+O6bzGkLVCSUB9oPsM7XZFFEQZUViIFFg7n0dFkFlomu9JgIfh9Nem+kEYQuWzYNw2hFiqf5xm6Ki1IWWBIF4RPBw8lPH4H//3sdaTs656GKayOHExOm7qiMSMRADc4lEK3w5LEuwWrNzHzyrmf7pUMClBuyT275d/15vkSBEBYxS52sl2eOtImICiMJlyOqfyKR+VAbmJCZOcjdi2MLdlQeGQaWwhKYke3wY7oQB9GwebRS5zQcoCQbwiCAQC5cJy+pl2+WZ1x8fysOzKMmQWZ2LaqWkAgL9v/q3sax7U3Cz3MISi7Hem2ISFvZSnBU/1+p1ki7ONn8zeRbsttXzRIOqhr4DtZfqzBdSjIRQ1HDQx1WnSEArHxnNJxv3N6bMs9G9S3ea+V1IWCOIVwoVn8RpQdwAAVUVIcyGRs1sAS64sMeu8xuDtxOZyMGRZ4KOxf2N8d/I7nf0mKQuBUYbHmIi6srCie9nzNFQmOkaYZpJ3dlApC6PaaltyFCGT5uTIHdMsc2INy0J1L1WZdVtzbgRIWSCIVwpnO2fO+c/tflYWXNqbtBc7H+ws1/w3Xt5QHluzbHJ5lIU8SZ7e/u9Pfo+55+caN1n9102+vyHUlQV3B/3VEasKzcJN22pR34a4mpoDf3fu36K5tyAAID4qyKTxz3O5ScM+7VbPnOKYHVIWCOIVoqFfQ855h5AOCHRR1cX45sQ3mpeYxEeHPlIeCwVCq6Vc9nFkFxeFsiBn5Lwlqxv6NtRqU68iWt+nPu/8a26uMU4QkR3w+u/ctviZ/GN1kJKXgnf3vIvDjw4D4DptKoqFEVzUlYWbT3OwbjQ3F4UllIVm4arMpA+e5xsc/+t/9znn6g6X6jlcbAVSFgjiFeLDmA85524ObryhjDdf3sTF9IvovKkzoldHc6Im5IxcZzEq9bnuZt1F879VPgp/dvuzvOIbjcKyUCwrxq6HuxCzJgZdNncxKvGUwjri5+yHxV0Ww9/ZH4MjB+Ojxh8ZuFIHTYdyz1uZNs/009Nx9flVTEyYCECVxnpc43E6s4u+6qhvQ3zYobZWuWe+2hHlRd1PIuGO/lTOfDjwFL+yJSymLMyYMQOtW7eGi4sLvLy8eMc8fvwYvXv3houLCwICAvDFF1+gpKSEdyxBEOXHx8kHV967gu9bfY/NfdgCRpoVN/MkeXh719sYvm84nhexP3rD9w1HjjgHAPDb5d/QbUs3rLy+Umt+fSb8JgFNzPU2DOJq76qM/FDkJHhR9ALvH3xfqTBkFGbg+svrAIDXar2mvFYRveHj5AN/F38cHngYk1tOxrAGw8ov2MgDbJ0IE3hR+EJ5nCfJU0ZqaG4pVTXKs22vbjlwtBfBXsOZ0dxhkwDgpKaAlBiZxVEdS8hkTiwmnUQiwcCBAzF27FjefplMht69e0MikeDUqVNYvXo1Vq1ahSlTplhKJIIgwG4PvBXxFiJ9IgEA7vbcfe9Om/gzV355jK0zsfvhbgDA/Ivz8fnRz5X9V59fRWo+T6hg6T2c7CqujLJAIIC90F6r/cyzM0jKTQIA/HLhF2X76OjR6BbWjTNWETmicIRzsXfBnv57lP0mlfoeuR944/+AGqalZhbLxJxcEPez7ysdLKu+slB2bUH9WgG0EyAlvSgo89y6UK+SqWnJ0ITPwuUgEmFkm5pwtBPio451zC5febGYsjB9+nR8+umniI7mT4164MAB3Lx5E2vXrkXjxo3Rs2dP/Pjjj1i8eDEkEv2JVAiCMB8CgQDb+21Xnos1qyOWcurpKQBcB7v9yftxMvUkdj7YiZOpJ3Xe4+DAgzr7LIWmf4aCIbuHAOBGf7jau2JKK+6DimIrQx11hafblm5aCZJ0UqMV0Mi0ypAl8hI0W9uM07bg4gLl56yw9FRVBjYLQWSQOz7pUrdc84R4Oxud+6A8qCsotfxd9Y7ly/JoLxJgSp8GuD49HuF++q+3Blaze5w+fRrR0dEIDFSZQOPj45Gbm4sbN27ovE4sFiM3N5fzIgiifNTyrGXUOIZhtDI9fnjoQ3xz4htOiOSQ+kOUx5NiJ8HVvuJ//ELcQ3jb86TsVom65UEAATwcPTjj+J7cNct4/3b5t/KKqZMtd7WLYF3OuKw8VmTFrKp4ONlj38T2ZY4SWDe6JWb2j0aTGt6oVUGLb9MaXgC0EzNJZXIsPnIfV1KyleeaKLYhNLdMbAWrSZWWlsZRFAAoz9PSdJv3Zs2aBU9PT+UrNDTUonISBKEiozBDmT9BF9Vcq2Fw5GDluXr65YpEPXOkJhKZRFkEDGDrZ2gmYtL15K6uWElkEk6RLHNy/cV1vf3qxcsIbVrX8cPgljUAAMIKKuGoKE5VoqEsbLqQgrn776DvYtYqxGfpqFI+C19//TUEAoHe1+3btw1PVA4mT56MnJwc5Sslxcg87QRBlJsdD3Yoj9sEt+Ed83WLrzn1MlzseLIYVgDqTothHmFYGa9yyPzr+l+csXwKja4cDV6OXsrjGy9vYMyBMeWUlB9D4XNNA5ta5L5E2VFkhdS0HBy7y42O4HOANHdGSXNjUtXJzz77DMOHD9c7plYt48yZQUFBOHfuHKctPT1d2acLR0dHODpaL9kLQbxK9KnVB+mF6TiXxv5f/fXyr8q+7+O+R49/eijPV/VYhcsZl9EhpAMYMBAKhJAzcvg6+1a43ACbQ0JBtF80mgU1g7uDO/IkeVicuFjZt/X1rbzX60p/rVmZ80L6BTNIq42++hQVUWOjKtOtQaDhQWVAUUxKXVnILpRg/w3u31JltCyYpCz4+/vD3988lbDi4uIwY8YMZGRkICCAjc0+ePAgPDw80KBBA7PcgyAI4wlxC8GT/CcAgC19tuBl0Us0DWwKoUCI2LWxnLG1PGuhult1TltsYCxiA1XjlnZdikJpIWL8YywvPA8CgQDre6/Hv/f/xcdNPwbAFtPSDO9Ufx+h7qHKqpJfNP+Cd151y4ICqUwKe5F29EVZUHjKvyx6qXNMVY+EsDSRQZbJfGmvtCyolIGJGxO1xvElbapSyoIpPH78GJmZmXj8+DFkMhkSExMBAHXq1IGbmxu6d++OBg0aYOjQoZgzZw7S0tLw3XffYdy4cWQ5IAgrMKbRGEw9NRWv134dET4RRl3zSdNPsOjSInzRTHthbR3c2twimkxDv4acqAgPBw+kghveqb7wLu++HDse7ED/Ov0R5Mpv4dS0LABAYUkhPEXa7aaiyHGhWQY70CWQY+kgZaF8WCo08dR9NifGvAN3lP4SfAmadl55qtVmq46NCiymLEyZMgWrV69WnjdpwiZkOXLkCDp27AiRSIRdu3Zh7NixiIuLg6urK4YNG4YffvjBUiIRBKGHN+q+gbpedRHmGWZwrACsuXVUw1F4rdZrOhdWW4OvloJ6yFt1t+oYG8OfG0YBn7KQLc7mbTeVbfe2aSkKgHadDfUaHITpqGd4NCcFEhkAILNAvxPw/hvaTvyWSEFtTiwm3apVq8AwjNarY8eOyjFhYWHYs2cPCgsL8fz5c8ybNw92dhbTXwiCMEC0fzQ8HDwMjlOkGRYIBJVGUQBg1HszBN9T/ZJE81TWLGH4Iys0y1ITptOjtNBTVHD5/wbKg1QmR1ahVKvdxcG2v2PbVmUIgrAJLg65yDmvrNUO1aM0gLKZ8/m2V/Yk7eEZaTp8xa7GRI/R8p+o612+REWvIrPfbITpr0dh1YgWVpXj0ctCzvm+ie1waFIHg1kfrQ0pCwRBGMRB5MCpTllZ98wTUhI45zv7mV6Su6ZnTfzd62/zCKQBXxrgYVHD0LZ6W/zz+j/Ktnnt51nk/lUZT2d7DGsdrlWu2pL8cfSBVpsiMZOCyCAP1Alw0xpna9i23YMgCJtBPWmRZvXKyoJ6yezR0aO1imgZSyP/RpjXYR4e5T7Cb5d/g4PQAQzDlKueAaCqKKmOwheinnc9/NjmR7jau6KWl3Eh6kTF8l5cGNacfqQ8n7VXO++Qt6sqamZyz8gKkcsckGWBIAijGB09GgCbjMla4ZDlZVKzScrjmp41yzVXfHi8Mq21RC7B50c/L3c2R0WlTAXreq3jnPer00+r4BVhOyj8IurqsRQcua2KjoiuXn6n2IqCLAsEQRjFwHoDEeMfY3QdCVsk2k9V2E5XBkpTUN+OOfDoAJrdbYZBkYPKPJ+XkxfnPNqfvxAfYZsociVIZHLeLSWAGylR3bvybOeRZYEgCKMQCASI8IkwW/Iha9DQryEWd1mM/QP2myWzpEAg4ChPux7uKtd8EplqIZncYnK55iIqHkc7NkpIUiLXWekyq5D9jqt7OSPM1/aqS+qClAWCIF4p2oe0R7BbsNnm2/jaRrQIYj3s/ZzKVjQrqzgLaQVpKJAWAAB61+qNwfUHG7iKsDXs7VTpniUa9R9q+LA1UsQlbHvLmj4VK1w5oW0IgiCIcuBk54S+dfriXNo5iGVik68vkBagxz89UFhSiAhvNnOmj1PlWkgIFkUWRkmJHGKpTNnu5+aIIE8nPM4sRH4x69dSXmfYioYsCwRBEOVEkWGxqKQIMrmMN1+CLu5l3UNhCRt7fyfrDgDAzd72Q+kIbRzUakMUSlTKwu+Dm+BcUiYA4E46W5vkn0tPKl7AckDKAkEQRDlRRDEUy4rx7p53MXDnQMjkMgNXsaQVaKf+zSjMMKt8RMWg7uC49ZKqBkmrWtapvGpOSFkgCIIoJ052rLKQXZyNGy9v4G7WXTzN1y4WxAdfKexwj3BzikdUEIptCJmcwaFb3O81Poqb06NLZECFyWUOSFkgCIIoJ4ptiIKSAmVbvlS7DDEfhdJCrba44DjzCEZUKOplpq+l5nD6NLM0xtWuXNYGUhYIgiDKiUJZyBGrFghjlQWJXLtCobElwgnbwl6k22mxQMzdlrL1ktSaVC5pCYIgbBC+zI25klyjrt10Z5O5xSGshL1Qe0n1dXUAAFx5ks0dS8oCQRDEq4WDyEGrTd3KoA9NpaJPrT5mkYmoeIQ8lSNflmZsdCpN2KRAnxXCFqE8CwRBEOWknnc9rTZTcy6MiR6DCJ8IdKnRxVxiETaEhzN3uVX3b6gMVC5pCYIgbBCBQIB21dtx2rLF2SbNEeQahPjweNgJ6RmuMuNsL+Jt93bhWp9oG4IgCOIVpLpbdc75ksQlOosJKVBP3vSy+KVF5CIqliIp15FRYUEYGhfGaSdlgSAI4hVEs2IkAGVmRl2oJ26q61XX3CIRNsCcAY0AAFHB3HLUlc1ngZQFgiAIM9C3dl+ttnyJ/vBJqVyqPG4d3NrsMhHWZdmQWPRrorI4jWpbU3nsQJYFgiCIV48Q9xB82fxLTpui1oMm6QXpOJ92Xml5EECgzAJJVB183bh+Co5qTo3k4EgQBPGKMqT+EGx4bYPyXJdlofs/3TFy/0icTD0JgE0XLRTQz3FV4M+hscpjO41QSnUFwctFO9zWlqG/ToIgCDMhEAgQ5RulPP/mxDe84xSOjUefHAUAONs5W144okJoFq4qLy7Soyy4OvJHTdgqpCwQBEFYCBmjXXny6vOryuNDjw4BADKLMytMJsKyqDsuCsBVFjJyVbk3XOwrV4gsKQsEQRAVBMMwGH94vOoc+kMricqHekikQCPgYXuiqmy1k0PlWn4tJm1ycjJGjRqFmjVrwtnZGbVr18bUqVMhkXCLply9ehXt2rWDk5MTQkNDMWfOHEuJRBAEUSHU9WbDIBUFphQcSTmCLHGWNUQiKghNPwV11HMwVLZoCIvZQW7fvg25XI4//vgDderUwfXr1zFmzBgUFBRg3rx5AIDc3Fx0794dXbt2xbJly3Dt2jWMHDkSXl5eeP/99y0lGkEQhEXpENIB97LuaaV8flbwjHd8iFtIRYhFVADqfgoeTvacPn93R6RkFgFg/VsqExZTFnr06IEePXooz2vVqoU7d+5g6dKlSmXh77//hkQiwV9//QUHBwdERUUhMTER8+fPJ2WBIIhKy7Enx5THEplEWWhKcw9bwZP8JxUiF2F5BAIB5g2MQXahBDV8XTh9Wz5sjZYzD6NxqJd1hCsHFWoHycnJgY+PylP09OnTaN++PRwcVCEk8fHxuHPnDrKy+E11YrEYubm5nBdBEIQtoZ5v4X72feUxX3VKourxZmwIRrerpdUe6OGE5J97499xbawgVfmoMGXh/v37+O233/DBBx8o29LS0hAYGMgZpzhPS0vjnWfWrFnw9PRUvkJDQy0nNEEQRBloWa2l8vjtXW/jcsZlAMDLIv76D/M7zq8QuQiirJisLHz99dcQCAR6X7dv3+Zck5qaih49emDgwIEYM2ZMuQSePHkycnJylK+UlJRyzUcQBGFpPjj4ARiGwe+Jv2v1hXuEo1tYNytIRRDGY7LPwmeffYbhw4frHVOrlsr88vTpU3Tq1AmtW7fGn3/+yRkXFBSE9PR0TpviPCgoiHduR0dHODo68vYRBEHYIkUlRZDIuZFgrYNbY3T0aDQLbGYlqQjCeExWFvz9/eHv72/U2NTUVHTq1AmxsbFYuXIlhEKuISMuLg7ffvstpFIp7O1Zr9GDBw8iIiIC3t7epopGEARhM4xsOBJ/Xf9LeV5cUszpr+NVB82Dmle0WARRJizms5CamoqOHTuiRo0amDdvHp4/f460tDSOL8LgwYPh4OCAUaNG4caNG9i4cSMWLVqESZMmWUosgiCICmFog6Gc890PdyuPu4V1w/uNKOKLqDwIGIaxSAqxVatWYcSIEbx96re8evUqxo0bh/Pnz8PPzw8TJkzAV199ZfR9cnNz4enpiZycHHh4eJRbboIgCHOx9MpSLElcotV+bdg1K0hDENoYu4ZaLM/C8OHDDfo2AECjRo1w/PhxS4lBEARhNTKLqOYDUTWoXPkmCYIgKhFOdk5abdVcq1lBEoIoH6QsEARBWAh7ob1Wm66UzwRhy5CyQBAEYSE8HMiPiqgakLJAEARhIRr5N9Jq83Hy4RlJELYNKQsEQRAWomlgU622RZ0WWUESgigfpCwQBEFYkEtDLnHOGwc0to4gBFEOSFkgCIKwIPYilZPj580+t6IkBFF2SFkgCIKoIKL9oq0tAkGUCYslZSIIgiBY/uj6Bx7lPeL1YSCIygApCwRBEBamdfXWaI3W1haDIMoMbUMQBEEQBKEXUhYIgiAIgtALKQsEQRAEQeiFlAWCIAiCIPRCygJBEARBEHohZYEgCIIgCL2QskAQBEEQhF5IWSAIgiAIQi+kLBAEQRAEoZdKn8GRYRgAQG5urpUlIQiCIIjKhWLtVKyluqj0ykJeXh4AIDQ01MqSEARBEETlJC8vD56enjr7BYwhdcLGkcvlePr0Kdzd3SEQCMwyZ25uLkJDQ5GSkgIPDw+zzGltqtp7qmrvB6D3VFmg91Q5oPdkHAzDIC8vD8HBwRAKdXsmVHrLglAoREhIiEXm9vDwqDJ/ZAqq2nuqau8HoPdUWaD3VDmg92QYfRYFBeTgSBAEQRCEXkhZIAiCIAhCL6Qs8ODo6IipU6fC0dHR2qKYjar2nqra+wHoPVUW6D1VDug9mZdK7+BIEARBEIRlIcsCQRAEQRB6IWWBIAiCIAi9kLJAEARBEIReSFkgCIIgCEIvr4SysHjxYoSHh8PJyQktW7bEuXPn9I7fvHkzIiMj4eTkhOjoaOzZs4fTzzAMpkyZgmrVqsHZ2Rldu3bFvXv3LPkWtDDlPS1fvhzt2rWDt7c3vL290bVrV63xw4cPh0Ag4Lx69Ohh6bfBwZT3tGrVKi15nZycOGMq2/fUsWNHrfckEAjQu3dv5Rhrfk/Hjh1Dnz59EBwcDIFAgH///dfgNQkJCWjatCkcHR1Rp04drFq1SmuMqf8/zYmp72nr1q3o1q0b/P394eHhgbi4OOzfv58zZtq0aVrfUWRkpAXfBRdT31NCQgLv311aWhpnXGX6nvj+nwgEAkRFRSnHWPt7mjVrFpo3bw53d3cEBASgX79+uHPnjsHrrLU+VXllYePGjZg0aRKmTp2KS5cuISYmBvHx8cjIyOAdf+rUKQwaNAijRo3C5cuX0a9fP/Tr1w/Xr19XjpkzZw5+/fVXLFu2DGfPnoWrqyvi4+NRXFxsk+8pISEBgwYNwpEjR3D69GmEhoaie/fuSE1N5Yzr0aMHnj17pnytX7++It4OANPfE8BmMVOX99GjR5z+yvY9bd26lfN+rl+/DpFIhIEDB3LGWet7KigoQExMDBYvXmzU+KSkJPTu3RudOnVCYmIiJk6ciNGjR3MW17J87+bE1Pd07NgxdOvWDXv27MHFixfRqVMn9OnTB5cvX+aMi4qK4nxHJ06csIT4vJj6nhTcuXOHI3NAQICyr7J9T4sWLeK8l5SUFPj4+Gj9X7Lm93T06FGMGzcOZ86cwcGDByGVStG9e3cUFBTovMaq6xNTxWnRogUzbtw45blMJmOCg4OZWbNm8Y5/6623mN69e3PaWrZsyXzwwQcMwzCMXC5ngoKCmLlz5yr7s7OzGUdHR2b9+vUWeAfamPqeNCkpKWHc3d2Z1atXK9uGDRvG9O3b19yiGo2p72nlypWMp6enzvmqwve0YMECxt3dncnPz1e2Wft7UgCA2bZtm94xX375JRMVFcVpe/vtt5n4+HjleXk/I3NizHvio0GDBsz06dOV51OnTmViYmLMJ1g5MOY9HTlyhAHAZGVl6RxT2b+nbdu2MQKBgElOTla22dL3xDAMk5GRwQBgjh49qnOMNdenKm1ZkEgkuHjxIrp27apsEwqF6Nq1K06fPs17zenTpznjASA+Pl45PikpCWlpaZwxnp6eaNmypc45zUlZ3pMmhYWFkEql8PHx4bQnJCQgICAAERERGDt2LF6+fGlW2XVR1veUn5+PsLAwhIaGom/fvrhx44ayryp8TytWrMA777wDV1dXTru1vidTMfR/yRyfkbWRy+XIy8vT+r907949BAcHo1atWnj33Xfx+PFjK0loPI0bN0a1atXQrVs3nDx5UtleFb6nFStWoGvXrggLC+O029L3lJOTAwBaf0vqWHN9qtLKwosXLyCTyRAYGMhpDwwM1NqPU5CWlqZ3vOJfU+Y0J2V5T5p89dVXCA4O5vxB9ejRA2vWrMHhw4cxe/ZsHD16FD179oRMJjOr/HyU5T1FRETgr7/+wvbt27F27VrI5XK0bt0aT548AVD5v6dz587h+vXrGD16NKfdmt+Tqej6v5Sbm4uioiKz/C1bm3nz5iE/Px9vvfWWsq1ly5ZYtWoV9u3bh6VLlyIpKQnt2rVDXl6eFSXVTbVq1bBs2TL8888/+OeffxAaGoqOHTvi0qVLAMzzm2NNnj59ir1792r9X7Kl70kul2PixIlo06YNGjZsqHOcNdenSl91kjCNn3/+GRs2bEBCQgLHIfCdd95RHkdHR6NRo0aoXbs2EhIS0KVLF2uIqpe4uDjExcUpz1u3bo369evjjz/+wI8//mhFyczDihUrEB0djRYtWnDaK9v3VJVZt24dpk+fju3bt3P293v27Kk8btSoEVq2bImwsDBs2rQJo0aNsoaoeomIiEBERITyvHXr1njw4AEWLFiA//3vf1aUzDysXr0aXl5e6NevH6fdlr6ncePG4fr16xXqM2EqVdqy4OfnB5FIhPT0dE57eno6goKCeK8JCgrSO17xrylzmpOyvCcF8+bNw88//4wDBw6gUaNGesfWqlULfn5+uH//frllNkR53pMCe3t7NGnSRClvZf6eCgoKsGHDBqN+sCryezIVXf+XPDw84OzsbJbv3Vps2LABo0ePxqZNm7TMwpp4eXmhXr16Nvkd6aJFixZKeSvz98QwDP766y8MHToUDg4Oesda63saP348du3ahSNHjiAkJETvWGuuT1VaWXBwcEBsbCwOHz6sbJPL5Th8+DDnqVSduLg4zngAOHjwoHJ8zZo1ERQUxBmTm5uLs2fP6pzTnJTlPQGsh+yPP/6Iffv2oVmzZgbv8+TJE7x8+RLVqlUzi9z6KOt7Ukcmk+HatWtKeSvr9/T/7d17XFVV3sfxLyoHTAVEESSF0EzNvKSmUjaaoWTmaPKU+phpmvZy0FKmi8wzDmJTWlN5adDKl6FZ6qSTWpmaN+wZAzMvk6YyangpBbvIRU00Wc8fPZyZI7ABZXM4+Hm/XvuVZ+11Fr/FAs6vtdfaW/p1a1R+fr4eeeSRUr9OZY5TeZX2u1QR4+4OS5cu1WOPPaalS5e6bGstydmzZ3XkyJEqOUYl2bNnjzNeTx0n6dcdB4cPHy5T4l3Z42SM0fjx47Vy5Upt3rxZERERpb7HrZ9P17Q80gMsW7bM+Pj4mIULF5r9+/ebsWPHmoCAAJOZmWmMMWb48OFm8uTJzvrbtm0ztWrVMq+88oo5cOCASUhIMN7e3mbv3r3OOjNmzDABAQFm9erV5quvvjIDBgwwERER5ueff66SfZoxY4ZxOBxmxYoV5tSpU84jLy/PGGNMXl6eefrpp01qaqrJyMgwGzduNB07djQtWrQwFy5cqJJ9SkxMNOvXrzdHjhwxO3fuNEOGDDG+vr7m66+/dum3J41Toe7du5vBgwcXKXf3OOXl5Zndu3eb3bt3G0nmtddeM7t37zbHjh0zxhgzefJkM3z4cGf9b775xtxwww3mmWeeMQcOHDBJSUmmZs2aZt26dc46pX2Pqlqf3nvvPVOrVi2TlJTk8ruUnZ3trPP73//epKSkmIyMDLNt2zYTFRVlGjZsaE6fPl0l+zRz5kyzatUqc+jQIbN3717z1FNPmRo1apiNGzc663jaOBV65JFHTNeuXYtt093jNG7cOOPv729SUlJcfpbOnz/vrFOVPp+qfbJgjDGvv/66CQsLMw6Hw3Tp0sWkpaU5z/Xo0cOMGDHCpf77779vbrnlFuNwOEybNm3MmjVrXM4XFBSYKVOmmODgYOPj42Puvfdek56eXhldcSpPn8LDw42kIkdCQoIxxpjz58+bPn36mKCgIOPt7W3Cw8PNmDFjKu0PwdX0aeLEic66wcHB5v777ze7du1yac/TxskYYw4ePGgkmU8//bRIW+4ep8ItdlcehX0YMWKE6dGjR5H3dOjQwTgcDtOsWTOTnJxcpF2r75HdytunHj16WNY35tftoY0bNzYOh8PceOONZvDgwebw4cNVtk8vvfSSad68ufH19TWBgYGmZ8+eZvPmzUXa9aRxMubXLYO1a9c2b731VrFtunuciuuPJJffkar0+cQjqgEAgKVqvWYBAABcO5IFAABgiWQBAABYIlkAAACWSBYAAIAlkgUAAGCJZAEAAFgiWQAAAJZIFgAAgCWSBQAAYIlkAQAAWCJZAAAAlkgWAFwTLy8vTZ061d1huFi8eLFatWolb29vBQQEuDscwOORLABV0MKFC+Xl5eU8fH19FRoaqujoaM2ZM0d5eXnuDrFEn3/+uaZOnars7Gy3fP2DBw9q5MiRat68uebPn6+33nqrxLqffPJJlUt0gKqolrsDAFCyadOmKSIiQpcuXVJmZqZSUlI0ceJEvfbaa/rwww/Vrl07d4eon3/+WbVq/ftPyeeff67ExESNHDnSLf9Xn5KSooKCAs2ePVs333yzZd1PPvlESUlJJAxAKUgWgCqsb9++6ty5s/N1fHy8Nm/erAceeEC//e1vdeDAAdWuXduNEUq+vr5u/fpXOn36tCRx+QGoQFyGADxMr169NGXKFB07dkzvvvuuy7mDBw/qv/7rvxQYGChfX1917txZH374oUudwksc27ZtU1xcnIKCglSnTh09+OCD+v77713qfvnll4qOjlbDhg1Vu3ZtRUREaNSoUS51/nPNwtSpU/XMM89IkiIiIpyXUY4ePaoePXqoffv2xfapZcuWio6OLrXvc+fOVZs2beTj46PQ0FDFxsa6XO646aablJCQIEkKCgqyXE8xcuRIJSUlOftQeBRatmyZOnXqpHr16snPz09t27bV7NmzS40RqI5IFgAPNHz4cEnSp59+6iz7+uuv1a1bNx04cECTJ0/Wq6++qjp16mjgwIFauXJlkTYmTJigf/7zn0pISNC4ceP00Ucfafz48c7zp0+fVp8+fXT06FFNnjxZr7/+uoYNG6a0tLQS4xo0aJCGDh0qSZo5c6YWL16sxYsXKygoSMOHD9dXX32lffv2ubxnx44d+te//qVHHnnEss9Tp05VbGysQkND9eqrryomJkZvvvmm+vTpo0uXLkmSZs2apQcffFCSNG/ePC1evFiDBg0qtr0nnnhCvXv3liRnnIsXL5YkbdiwQUOHDlX9+vX10ksvacaMGerZs6e2bdtmGSNQbRkAVU5ycrKRZHbs2FFiHX9/f3P77bc7X997772mbdu25sKFC86ygoICc+edd5oWLVoUaTsqKsoUFBQ4yydNmmRq1qxpsrOzjTHGrFy5stQYjDFGkklISHC+/stf/mIkmYyMDJd62dnZxtfX1zz33HMu5U8++aSpU6eOOXv2bIlf4/Tp08bhcJg+ffqYy5cvO8v/+te/Gknm7bffdpYlJCQYSeb777+3jNsYY2JjY01xfwafeuop4+fnZ3755ZdS2wCuB8wsAB6qbt26zl0RP/30kzZv3qyHH35YeXl5+uGHH/TDDz/oxx9/VHR0tA4dOqTvvvvO5f1jx451mXa/++67dfnyZR07dkzSv6/5f/zxx87/c78W/v7+GjBggJYuXSpjjCTp8uXL+tvf/qaBAweqTp06Jb5348aNunjxoiZOnKgaNf79Z2vMmDHy8/PTmjVrrjm+/xQQEKBz585pw4YNFdou4KlIFgAPdfbsWdWrV0+SdPjwYRljNGXKFAUFBbkchdfwCxf+FQoLC3N5Xb9+fUnSmTNnJEk9evRQTEyMEhMT1bBhQw0YMEDJycnKz8+/6pgfffRRHT9+XP/7v/8r6dckICsry3lZpSSFCUzLli1dyh0Oh5o1a+Y8X1F+97vf6ZZbblHfvn3VpEkTjRo1SuvWravQrwF4EnZDAB7o22+/VU5OjnNrYEFBgSTp6aefLnGh4JXbCGvWrFlsvcL/6/fy8tKKFSuUlpamjz76SOvXr9eoUaP06quvKi0tTXXr1i133NHR0QoODta7776r3/zmN3r33XcVEhKiqKiocrdlp0aNGmnPnj1av3691q5dq7Vr1yo5OVmPPvqoFi1a5O7wgErHzALggQoX4hUmBs2aNZMkeXt7KyoqqtijcBaivLp166YXXnhBX375pd577z19/fXXWrZsWYn1//PSxpVq1qyp//7v/9aKFSt05swZrVq1SkOHDi0xcSkUHh4uSUpPT3cpv3jxojIyMpzny8sqVofDof79+2vu3Lk6cuSInnjiCb3zzjs6fPjwVX0twJORLAAeZvPmzXr++ecVERGhYcOGSfr1/4R79uypN998U6dOnSryniu3RJbFmTNnnLMMhTp06CBJlpciCtcelHQHx+HDh+vMmTN64okndPbs2VJ3QUhSVFSUHA6H5syZ4xLTggULlJOTo379+pXaRnli/fHHH11e16hRw3kDrGu5DAN4Ki5DAFXY2rVrdfDgQf3yyy/KysrS5s2btWHDBoWHh+vDDz90uSFSUlKSunfvrrZt22rMmDFq1qyZsrKylJqaqm+//Vb//Oc/y/W1Fy1apLlz5+rBBx9U8+bNlZeXp/nz58vPz0/3339/ie/r1KmTJOl//ud/NGTIEHl7e6t///7OD+bbb79dt912m5YvX67WrVurY8eOpcYSFBSk+Ph4JSYm6r777tNvf/tbpaena+7cubrjjjvKlHBYxfrkk08qOjpaNWvW1JAhQ/T444/rp59+Uq9evdSkSRMdO3ZMr7/+ujp06KDWrVtf1dcCPJpb92IAKFbh9sbCw+FwmJCQENO7d28ze/Zsk5ubW+z7jhw5Yh599FETEhJivL29zY033mgeeOABs2LFiiJtX7klcsuWLUaS2bJlizHGmF27dpmhQ4easLAw4+PjYxo1amQeeOAB8+WXX7q8T1dsnTTGmOeff97ceOONpkaNGsVuo3z55ZeNJPPiiy+W6/vy17/+1bRq1cp4e3ub4OBgM27cOHPmzBmXOuXZOvnLL7+YCRMmmKCgIOPl5eXcRrlixQrTp08f06hRI+NwOExYWJh54oknzKlTp8oVL1BdeBlzxTwjANhs9uzZmjRpko4ePVpkVwaAqodkAUClMsaoffv2atCggbZs2eLucACUAWsWAFSKc+fO6cMPP9SWLVu0d+9erV692t0hASgjZhYAVIqjR48qIiJCAQEB+t3vfqcXXnjB3SEBKCOSBQAAYIn7LAAAAEskCwAAwJLHL3AsKCjQyZMnVa9ePctbtwIAAFfGGOXl5Sk0NNTlia5X8vhk4eTJk2ratKm7wwAAwGOdOHFCTZo0KfG8xycLhQ/HOXHihPz8/NwcDQAAniM3N1dNmzYt9UFzHp8sFF568PPzI1kAAOAqlHYZnwWOAADAEskCAACwRLIAAAAskSwAAABLJAsAAMASyQIAALBEsgAAACyRLAAAAEskCwAAwBLJAgAAsESyAAAALHn8syEAAFXTTZPXFCk7OqOfGyLBtWJmAQAAWCJZAAAAlrgMAQCAG1x5maYqX6JhZgEAAFgiWQAAAJZIFgAAgCWSBQAAYIlkAQAAWCJZAAAAlmxNFm666SZ5eXkVOWJjYyVJFy5cUGxsrBo0aKC6desqJiZGWVlZdoYEAECVdNPkNUWOqsLWZGHHjh06deqU89iwYYMk6aGHHpIkTZo0SR999JGWL1+urVu36uTJkxo0aJCdIQEAgHKy9aZMQUFBLq9nzJih5s2bq0ePHsrJydGCBQu0ZMkS9erVS5KUnJys1q1bKy0tTd26dbMzNAAAUEaVtmbh4sWLevfddzVq1Ch5eXlp586dunTpkqKiopx1WrVqpbCwMKWmppbYTn5+vnJzc10OAABgn0q73fOqVauUnZ2tkSNHSpIyMzPlcDgUEBDgUi84OFiZmZkltjN9+nQlJibaGCkAwJOVdq2/Kt9WuaqqtJmFBQsWqG/fvgoNDb2mduLj45WTk+M8Tpw4UUERAgCA4lTKzMKxY8e0ceNGffDBB86ykJAQXbx4UdnZ2S6zC1lZWQoJCSmxLR8fH/n4+NgZLgAA+A+VMrOQnJysRo0aqV+/f0/9dOrUSd7e3tq0aZOzLD09XcePH1dkZGRlhAUAAMrA9pmFgoICJScna8SIEapV699fzt/fX6NHj1ZcXJwCAwPl5+enCRMmKDIykp0QAACP50mPoC6N7cnCxo0bdfz4cY0aNarIuZkzZ6pGjRqKiYlRfn6+oqOjNXfuXLtDAgAA5WB7stCnTx8ZY4o95+vrq6SkJCUlJdkdBgAAuEqVtnUSAIDyqk5T+Z6MB0kBAABLJAsAAMASyQIAALBEsgAAACyRLAAAAEskCwAAwBJbJwEAVUZpT4yEezCzAAAALJEsAAAASyQLAADAEmsWAACoBJ68HoOZBQAAYIlkAQAAWCJZAAAAllizAADwGMVd9+ex1fZjZgEAAFgiWQAAAJa4DAEAuK5wKaP8mFkAAACWSBYAAIAlkgUAAGCJNQsAgDK58lp/VbnO78m3UfYUzCwAAABLJAsAAMCS7ZchvvvuOz333HNau3atzp8/r5tvvlnJycnq3LmzJMkYo4SEBM2fP1/Z2dm66667NG/ePLVo0cLu0AAAJWBqv2qoKpd+bJ1ZOHPmjO666y55e3tr7dq12r9/v1599VXVr1/fWefll1/WnDlz9MYbb2j79u2qU6eOoqOjdeHCBTtDAwAAZWTrzMJLL72kpk2bKjk52VkWERHh/LcxRrNmzdIf//hHDRgwQJL0zjvvKDg4WKtWrdKQIUPsDA8AAJSBrTMLH374oTp37qyHHnpIjRo10u2336758+c7z2dkZCgzM1NRUVHOMn9/f3Xt2lWpqanFtpmfn6/c3FyXAwAA2MfWmYVvvvlG8+bNU1xcnP7whz9ox44devLJJ+VwODRixAhlZmZKkoKDg13eFxwc7Dx3penTpysxMdHOsAEAKJfqvsbD1pmFgoICdezYUS+++KJuv/12jR07VmPGjNEbb7xx1W3Gx8crJyfHeZw4caICIwYAAFeyNVlo3Lixbr31Vpey1q1b6/jx45KkkJAQSVJWVpZLnaysLOe5K/n4+MjPz8/lAAAA9rE1WbjrrruUnp7uUvavf/1L4eHhkn5d7BgSEqJNmzY5z+fm5mr79u2KjIy0MzQAAFBGtq5ZmDRpku688069+OKLevjhh/XFF1/orbfe0ltvvSVJ8vLy0sSJE/XnP/9ZLVq0UEREhKZMmaLQ0FANHDjQztAAAEAZ2Zos3HHHHVq5cqXi4+M1bdo0RUREaNasWRo2bJizzrPPPqtz585p7Nixys7OVvfu3bVu3Tr5+vraGRoAACgj2+/g+MADD+iBBx4o8byXl5emTZumadOm2R0KAAC4Cjx1EgBwVarKrYhhPx4kBQAALJEsAAAASyQLAADAEmsWAOA6xHoDlAczCwAAwBLJAgAAsMRlCADwcFxSuHZ8D60xswAAACyRLAAAAEskCwAAwBJrFgCgiuH6edVzvY8JMwsAAMASyQIAALBEsgAAACyxZgEAgCtcuUbhesfMAgAAsESyAAAALHEZAgBQIZi6r76YWQAAAJZIFgAAgCWSBQAAYIlkAQAAWCJZAAAAlkgWAACAJVuThalTp8rLy8vlaNWqlfP8hQsXFBsbqwYNGqhu3bqKiYlRVlaWnSEBANzopslrXA54BttnFtq0aaNTp045j3/84x/Oc5MmTdJHH32k5cuXa+vWrTp58qQGDRpkd0gAAKAcbL8pU61atRQSElKkPCcnRwsWLNCSJUvUq1cvSVJycrJat26ttLQ0devWze7QAABAGdg+s3Do0CGFhoaqWbNmGjZsmI4fPy5J2rlzpy5duqSoqChn3VatWiksLEypqakltpefn6/c3FyXAwAA2MfWmYWuXbtq4cKFatmypU6dOqXExETdfffd2rdvnzIzM+VwOBQQEODynuDgYGVmZpbY5vTp05WYmGhn2ACASsK6Bc9ga7LQt29f57/btWunrl27Kjw8XO+//75q1659VW3Gx8crLi7O+To3N1dNmza95lgBAEDxKnXrZEBAgG655RYdPnxYISEhunjxorKzs13qZGVlFbvGoZCPj4/8/PxcDgAAYJ9KTRbOnj2rI0eOqHHjxurUqZO8vb21adMm5/n09HQdP35ckZGRlRkWAACwYOtliKefflr9+/dXeHi4Tp48qYSEBNWsWVNDhw6Vv7+/Ro8erbi4OAUGBsrPz08TJkxQZGQkOyEA4BqwDgAVzdZk4dtvv9XQoUP1448/KigoSN27d1daWpqCgoIkSTNnzlSNGjUUExOj/Px8RUdHa+7cuXaGBAAAysnWZGHZsmWW5319fZWUlKSkpCQ7wwAAANfA9psyAQD+7cpLBEdn9HNTJEDZ8SApAABgiWQBAABYIlkAAACWWLMAADZiGyOqA2YWAACAJZIFAABgiWQBAABYYs0CAADldL2tRWFmAQAAWCJZAAAAlrgMAQBudL1NZ8MzMbMAAAAskSwAAABLJAsAAMASyQIAALBEsgAAACyRLAAAAEskCwAAwBLJAgAAsESyAAAALJEsAAAASyQLAADAEskCAACwRLIAAAAsVVqyMGPGDHl5eWnixInOsgsXLig2NlYNGjRQ3bp1FRMTo6ysrMoKCQAAlEGlJAs7duzQm2++qXbt2rmUT5o0SR999JGWL1+urVu36uTJkxo0aFBlhAQAAMrI9mTh7NmzGjZsmObPn6/69es7y3NycrRgwQK99tpr6tWrlzp16qTk5GR9/vnnSktLszssAABQRrYnC7GxserXr5+ioqJcynfu3KlLly65lLdq1UphYWFKTU0tsb38/Hzl5ua6HAAAwD617Gx82bJl2rVrl3bs2FHkXGZmphwOhwICAlzKg4ODlZmZWWKb06dPV2JiYkWHCgAV4qbJa9wdAlDhbJtZOHHihJ566im999578vX1rbB24+PjlZOT4zxOnDhRYW0DAICibEsWdu7cqdOnT6tjx46qVauWatWqpa1bt2rOnDmqVauWgoODdfHiRWVnZ7u8LysrSyEhISW26+PjIz8/P5cDAADYx7bLEPfee6/27t3rUvbYY4+pVatWeu6559S0aVN5e3tr06ZNiomJkSSlp6fr+PHjioyMtCssAABQTrYlC/Xq1dNtt93mUlanTh01aNDAWT569GjFxcUpMDBQfn5+mjBhgiIjI9WtWze7wgIAAOVk6wLH0sycOVM1atRQTEyM8vPzFR0drblz57ozJAAAcIVKTRZSUlJcXvv6+iopKUlJSUmVGQYAACgHt84sAACqBrZ8wgoPkgIAAJZIFgAAgCWSBQAAYIk1CwBQxV25nuDojH5uigTXK2YWAACAJZIFAABgiWQBAABYIlkAAACWSBYAAIAlkgUAAGCJZAEAAFgiWQAAAJZIFgAAgCWSBQAAYInbPQNwKu4xxdxauOrhcdKobMwsAAAASyQLAADAEskCAACwRLIAAAAskSwAAABLJAsAAMASWycBWLpymx5bKYHrDzMLAADAEskCAACwZGuyMG/ePLVr105+fn7y8/NTZGSk1q5d6zx/4cIFxcbGqkGDBqpbt65iYmKUlZVlZ0gAAKCcbF2z0KRJE82YMUMtWrSQMUaLFi3SgAEDtHv3brVp00aTJk3SmjVrtHz5cvn7+2v8+PEaNGiQtm3bZmdYACCp9PUYrNcAfmVrstC/f3+X1y+88ILmzZuntLQ0NWnSRAsWLNCSJUvUq1cvSVJycrJat26ttLQ0devWzc7QAABAGVXamoXLly9r2bJlOnfunCIjI7Vz505dunRJUVFRzjqtWrVSWFiYUlNTS2wnPz9fubm5LgcAALCP7Vsn9+7dq8jISF24cEF169bVypUrdeutt2rPnj1yOBwKCAhwqR8cHKzMzMwS25s+fboSExNtjhqAp+MSAlBxbJ9ZaNmypfbs2aPt27dr3LhxGjFihPbv33/V7cXHxysnJ8d5nDhxogKjBQAAV7J9ZsHhcOjmm2+WJHXq1Ek7duzQ7NmzNXjwYF28eFHZ2dkuswtZWVkKCQkpsT0fHx/5+PjYHTYAAPh/lX6fhYKCAuXn56tTp07y9vbWpk2bnOfS09N1/PhxRUZGVnZYAACgBLbOLMTHx6tv374KCwtTXl6elixZopSUFK1fv17+/v4aPXq04uLiFBgYKD8/P02YMEGRkZHshABQ4a5cwwCg7GxNFk6fPq1HH31Up06dkr+/v9q1a6f169erd+/ekqSZM2eqRo0aiomJUX5+vqKjozV37lw7QwIAAOVka7KwYMECy/O+vr5KSkpSUlKSnWEAAIBrwFMnAaCMuJSB6xUPkgIAAJZIFgAAgCWSBQAAYIk1CwAqFLdZBqofZhYAAIAlkgUAAGCJZAEAAFhizQKAaoF7IAD2YWYBAABYIlkAAACWuAwBXKE6b/2riL5dT98fAL9iZgEAAFgiWQAAAJZIFgAAgCWSBQAAYIlkAQAAWCJZAAAAlkgWAACAJZIFAABgiWQBAABYIlkAAACWuN0zUEVV59sqA/AszCwAAABLJAsAAMCSrcnC9OnTdccdd6hevXpq1KiRBg4cqPT0dJc6Fy5cUGxsrBo0aKC6desqJiZGWVlZdoYFAADKwdY1C1u3blVsbKzuuOMO/fLLL/rDH/6gPn36aP/+/apTp44kadKkSVqzZo2WL18uf39/jR8/XoMGDdK2bdvsDA2ABynu0dGs4QAqj63Jwrp161xeL1y4UI0aNdLOnTv1m9/8Rjk5OVqwYIGWLFmiXr16SZKSk5PVunVrpaWlqVu3bnaGBwAAyqBS1yzk5ORIkgIDAyVJO3fu1KVLlxQVFeWs06pVK4WFhSk1NbXYNvLz85Wbm+tyAAAA+1Ta1smCggJNnDhRd911l2677TZJUmZmphwOhwICAlzqBgcHKzMzs9h2pk+frsTERLvDBZzYwnhtuIQAeL5Km1mIjY3Vvn37tGzZsmtqJz4+Xjk5Oc7jxIkTFRQhAAAoTqXMLIwfP14ff/yxPvvsMzVp0sRZHhISoosXLyo7O9tldiErK0shISHFtuXj4yMfHx+7QwYAAP/P1pkFY4zGjx+vlStXavPmzYqIiHA536lTJ3l7e2vTpk3OsvT0dB0/flyRkZF2hgYAAMrI1pmF2NhYLVmyRKtXr1a9evWc6xD8/f1Vu3Zt+fv7a/To0YqLi1NgYKD8/Pw0YcIERUZGshMCAIAqwtZkYd68eZKknj17upQnJydr5MiRkqSZM2eqRo0aiomJUX5+vqKjozV37lw7wwIAAOVga7JgjCm1jq+vr5KSkpSUlGRnKAAA4Crx1ElUK5WxzZGtgACuNzxICgAAWCJZAAAAlkgWAACAJdYsoEy45XH52PH9qow2K0tpX5efL6BqYWYBAABYIlkAAACWSBYAAIAl1iygyirL/Qzcdc39SlUlDnewo+/X8/cTqIqYWQAAAJZIFgAAgCWSBQAAYIlkAQAAWCJZAAAAlkgWAACAJZIFAABgiWQBAABYIlkAAACWuIMjPAp39gOAysfMAgAAsESyAAAALJEsAAAAS6xZQLEqY23AlV/jyidKAgCqBmYWAACAJZIFAABgydZk4bPPPlP//v0VGhoqLy8vrVq1yuW8MUZ/+tOf1LhxY9WuXVtRUVE6dOiQnSEBAIBysnXNwrlz59S+fXuNGjVKgwYNKnL+5Zdf1pw5c7Ro0SJFRERoypQpio6O1v79++Xr62tnaEClKsv6DDvWiXBfCgAVwdZkoW/fvurbt2+x54wxmjVrlv74xz9qwIABkqR33nlHwcHBWrVqlYYMGWJnaAAAoIzctmYhIyNDmZmZioqKcpb5+/ura9euSk1NLfF9+fn5ys3NdTkAAIB93LZ1MjMzU5IUHBzsUh4cHOw8V5zp06crMTHR1thQutKm1Yub/mZrJCoSl1iAyuNxuyHi4+OVk5PjPE6cOOHukAAAqNbcliyEhIRIkrKyslzKs7KynOeK4+PjIz8/P5cDAADYx23JQkREhEJCQrRp0yZnWW5urrZv367IyEh3hQUAAK5g65qFs2fP6vDhw87XGRkZ2rNnjwIDAxUWFqaJEyfqz3/+s1q0aOHcOhkaGqqBAwfaGRbgwlOvfXO7bACVxdZk4csvv9Q999zjfB0XFydJGjFihBYuXKhnn31W586d09ixY5Wdna3u3btr3bp13GMBAIAqxNZkoWfPnjLGlHjey8tL06ZN07Rp0+wMAwAAXAOeOolK447pfk+9xHA1rqe+AqhcHrd1EgAAVC6SBQAAYIlkAQAAWGLNQhXnKdvjuF4OANUXMwsAAMASyQIAALBEsgAAACyxZqEEdqwVcNf6A09Z9wAAqJqYWQAAAJZIFgAAgCUuQ1yDa53eL267YWltlGWLIpcZqj62mgLwJMwsAAAASyQLAADAEskCAACwxJqFKqYirmWXt42qcv28qsQBAHDFzAIAALBEsgAAACxxGaICeco0uqfECQCoGphZAAAAlkgWAACAJZIFAABgiTULZcR1fgDA9YqZBQAAYIlkAQAAWKoSyUJSUpJuuukm+fr6qmvXrvriiy/cHRIAAPh/bk8W/va3vykuLk4JCQnatWuX2rdvr+joaJ0+fdrdoQEAAFWBZOG1117TmDFj9Nhjj+nWW2/VG2+8oRtuuEFvv/22u0MDAABy826IixcvaufOnYqPj3eW1ahRQ1FRUUpNTS32Pfn5+crPz3e+zsnJkSTl5uZWaGwF+ecrtD0AAK5VRX/WFbZnjLGs59Zk4YcfftDly5cVHBzsUh4cHKyDBw8W+57p06crMTGxSHnTpk1tiREAgKrCf5Y97ebl5cnf37/E8x53n4X4+HjFxcU5XxcUFOinn35SgwYN5OXlVSFfIzc3V02bNtWJEyfk5+dXIW26W3XrU3Xrj0SfPAV98gz0qWyMMcrLy1NoaKhlPbcmCw0bNlTNmjWVlZXlUp6VlaWQkJBi3+Pj4yMfHx+XsoCAAFvi8/PzqzY/ZIWqW5+qW38k+uQp6JNnoE+ls5pRKOTWBY4Oh0OdOnXSpk2bnGUFBQXatGmTIiMj3RgZAAAo5PbLEHFxcRoxYoQ6d+6sLl26aNasWTp37pwee+wxd4cGAABUBZKFwYMH6/vvv9ef/vQnZWZmqkOHDlq3bl2RRY+VycfHRwkJCUUud3iy6tan6tYfiT55CvrkGehTxfIype2XAAAA1zW335QJAABUbSQLAADAEskCAACwRLIAAAAsXRfJQnkfgb18+XK1atVKvr6+atu2rT755BOX88YY/elPf1Ljxo1Vu3ZtRUVF6dChQ3Z2oYjy9Gn+/Pm6++67Vb9+fdWvX19RUVFF6o8cOVJeXl4ux3333Wd3N1yUp08LFy4sEq+vr69LHU8bp549exbpk5eXl/r16+es485x+uyzz9S/f3+FhobKy8tLq1atKvU9KSkp6tixo3x8fHTzzTdr4cKFReq48xH15e3TBx98oN69eysoKEh+fn6KjIzU+vXrXepMnTq1yBi1atXKxl64Km+fUlJSiv25y8zMdKnnSeNU3O+Jl5eX2rRp46zj7nGaPn267rjjDtWrV0+NGjXSwIEDlZ6eXur73PX5VO2ThfI+Avvzzz/X0KFDNXr0aO3evVsDBw7UwIEDtW/fPmedl19+WXPmzNEbb7yh7du3q06dOoqOjtaFCxeqZJ9SUlI0dOhQbdmyRampqWratKn69Omj7777zqXefffdp1OnTjmPpUuXVkZ3JF3do8r9/Pxc4j127JjLeU8bpw8++MClP/v27VPNmjX10EMPudRz1zidO3dO7du3V1JSUpnqZ2RkqF+/frrnnnu0Z88eTZw4UY8//rjLh6u7H1Ff3j599tln6t27tz755BPt3LlT99xzj/r376/du3e71GvTpo3LGP3jH/+wI/xilbdPhdLT011ibtSokfOcp43T7NmzXfpy4sQJBQYGFvldcuc4bd26VbGxsUpLS9OGDRt06dIl9enTR+fOnSvxPW79fDLVXJcuXUxsbKzz9eXLl01oaKiZPn16sfUffvhh069fP5eyrl27mieeeMIYY0xBQYEJCQkxf/nLX5zns7OzjY+Pj1m6dKkNPSiqvH260i+//GLq1atnFi1a5CwbMWKEGTBgQEWHWmbl7VNycrLx9/cvsb3qME4zZ8409erVM2fPnnWWuXucCkkyK1eutKzz7LPPmjZt2riUDR482ERHRztfX+v3qCKVpU/FufXWW01iYqLzdUJCgmnfvn3FBXYNytKnLVu2GEnmzJkzJdbx9HFauXKl8fLyMkePHnWWVaVxMsaY06dPG0lm69atJdZx5+dTtZ5ZKHwEdlRUlLOstEdgp6amutSXpOjoaGf9jIwMZWZmutTx9/dX165dS2yzIl1Nn650/vx5Xbp0SYGBgS7lKSkpatSokVq2bKlx48bpxx9/rNDYS3K1fTp79qzCw8PVtGlTDRgwQF9//bXzXHUYpwULFmjIkCGqU6eOS7m7xqm8SvtdqojvkbsVFBQoLy+vyO/SoUOHFBoaqmbNmmnYsGE6fvy4myIsuw4dOqhx48bq3bu3tm3b5iyvDuO0YMECRUVFKTw83KW8Ko1TTk6OJBX5WfpP7vx8qtbJgtUjsK+8HlcoMzPTsn7hf8vTZkW6mj5d6bnnnlNoaKjLD9R9992nd955R5s2bdJLL72krVu3qm/fvrp8+XKFxl+cq+lTy5Yt9fbbb2v16tV69913VVBQoDvvvFPffvutJM8fpy+++EL79u3T448/7lLuznEqr5J+l3Jzc/Xzzz9XyM+yu73yyis6e/asHn74YWdZ165dtXDhQq1bt07z5s1TRkaG7r77buXl5bkx0pI1btxYb7zxhv7+97/r73//u5o2baqePXtq165dkirmb447nTx5UmvXri3yu1SVxqmgoEATJ07UXXfdpdtuu63Eeu78fHL77Z5RuWbMmKFly5YpJSXFZUHgkCFDnP9u27at2rVrp+bNmyslJUX33nuvO0K1FBkZ6fKwsTvvvFOtW7fWm2++qeeff96NkVWMBQsWqG3bturSpYtLuaeNU3W2ZMkSJSYmavXq1S7X9/v27ev8d7t27dS1a1eFh4fr/fff1+jRo90RqqWWLVuqZcuWztd33nmnjhw5opkzZ2rx4sVujKxiLFq0SAEBARo4cKBLeVUap9jYWO3bt69S10yUV7WeWbiaR2CHhIRY1i/8b3narEhX06dCr7zyimbMmKFPP/1U7dq1s6zbrFkzNWzYUIcPH77mmEtzLX0q5O3trdtvv90ZryeP07lz57Rs2bIy/cGqzHEqr5J+l/z8/FS7du0KGXd3WbZsmR5//HG9//77RaaFrxQQEKBbbrmlSo5RSbp06eKM15PHyRijt99+W8OHD5fD4bCs665xGj9+vD7++GNt2bJFTZo0sazrzs+nap0sXM0jsCMjI13qS9KGDRuc9SMiIhQSEuJSJzc3V9u3b6+Ux2pf7WO9X375ZT3//PNat26dOnfuXOrX+fbbb/Xjjz+qcePGFRK3lYp4VPnly5e1d+9eZ7yeOk7Sr1uj8vPz9cgjj5T6dSpznMqrtN8lT31E/dKlS/XYY49p6dKlLttaS3L27FkdOXKkSo5RSfbs2eOM11PHSfp1x8Hhw4fLlHhX9jgZYzR+/HitXLlSmzdvVkRERKnvcevn0zUtj/QAy5YtMz4+PmbhwoVm//79ZuzYsSYgIMBkZmYaY4wZPny4mTx5srP+tm3bTK1atcwrr7xiDhw4YBISEoy3t7fZu3evs86MGTNMQECAWb16tfnqq6/MgAEDTEREhPn555+rZJ9mzJhhHA6HWbFihTl16pTzyMvLM8YYk5eXZ55++mmTmppqMjIyzMaNG03Hjh1NixYtzIULF6pknxITE8369evNkSNHzM6dO82QIUOMr6+v+frrr1367UnjVKh79+5m8ODBRcrdPU55eXlm9+7dZvfu3UaSee2118zu3bvNsWPHjDHGTJ482QwfPtxZ/5tvvjE33HCDeeaZZ8yBAwdMUlKSqVmzplm3bp2zTmnfo6rWp/fee8/UqlXLJCUlufwuZWdnO+v8/ve/NykpKSYjI8Ns27bNREVFmYYNG5rTp09XyT7NnDnTrFq1yhw6dMjs3bvXPPXUU6ZGjRpm48aNzjqeNk6FHnnkEdO1a9di23T3OI0bN874+/ublJQUl5+l8+fPO+tUpc+nap8sGGPM66+/bsLCwozD4TBdunQxaWlpznM9evQwI0aMcKn//vvvm1tuucU4HA7Tpk0bs2bNGpfzBQUFZsqUKSY4ONj4+PiYe++916Snp1dGV5zK06fw8HAjqciRkJBgjDHm/Pnzpk+fPiYoKMh4e3ub8PBwM2bMmEr7Q3A1fZo4caKzbnBwsLn//vvNrl27XNrztHEyxpiDBw8aSebTTz8t0pa7x6lwi92VR2EfRowYYXr06FHkPR06dDAOh8M0a9bMJCcnF2nX6ntkt/L2qUePHpb1jfl1e2jjxo2Nw+EwN954oxk8eLA5fPhwle3TSy+9ZJo3b258fX1NYGCg6dmzp9m8eXORdj1pnIz5dctg7dq1zVtvvVVsm+4ep+L6I8nld6QqfT7xiGoAAGCpWq9ZAAAA145kAQAAWCJZAAAAlkgWAACAJZIFAABgiWQBAABYIlkAAACWSBYAAIAlkgUAAGCJZAEAAFgiWQAAAJZIFgAAgKX/A5ubpBdsQEyyAAAAAElFTkSuQmCC", + "text/plain": [ + "
" + ] + }, + "metadata": {}, + "output_type": "display_data" + } + ], + "source": [ + "controller = diffrax.PIDController(\n", + " rtol=0,\n", + " atol=0.005,\n", + " pcoeff=0.2,\n", + " icoeff=0.5,\n", + " dcoeff=0,\n", + " dtmin=2**-12,\n", + " dtmax=0.25,\n", + ")\n", + "\n", + "solver = diffrax.SPaRK()\n", + "# solver = diffrax.HalfSolver(diffrax.GeneralShARK())\n", + "\n", + "sol_pid_spark = diffrax.diffeqsolve(\n", + " terms,\n", + " solver,\n", + " t0,\n", + " t1,\n", + " dt0,\n", + " y0,\n", + " args,\n", + " saveat=diffrax.SaveAt(steps=True),\n", + " stepsize_controller=controller,\n", + " max_steps=2**16,\n", + ")\n", + "accepted_steps = sol_pid_spark.stats[\"num_accepted_steps\"]\n", + "rejected_steps = sol_pid_spark.stats[\"num_rejected_steps\"]\n", + "print(\n", + " f\"Accepted steps: {accepted_steps}, Rejected steps: {rejected_steps},\"\n", + " f\" total steps: {accepted_steps + rejected_steps}\"\n", + ")\n", + "\n", + "# Plot the solution on ax1 and the density of ts on ax2\n", + "fig, (ax1, ax2) = plt.subplots(2, 1, figsize=(6, 8))\n", + "ax1.plot(sol_pid_spark.ts, sol_pid_spark.ys[:, 0], label=\"y_1\")\n", + "ax1.plot(sol_pid_spark.ts, sol_pid_spark.ys[:, 1], label=\"y_2\")\n", + "ax1.plot(sol_pid_spark.ts, sol_pid_spark.ys[:, 2], label=\"y_3\")\n", + "ax1.set_title(\"SDE solution\")\n", + "ax1.legend()\n", + "\n", + "# Plot the density of ts\n", + "# sol_pid.ts is padded with inf values at the end, so we remove them\n", + "padding_idx = jnp.argmax(jnp.isinf(sol_pid_spark.ts))\n", + "ts = sol_pid_spark.ts[:padding_idx]\n", + "ax2.hist(ts, bins=100)\n", + "ax2.set_title(\"Density of ts\")\n", + "\n", + "plt.show()" + ] + }, + { + "cell_type": "markdown", + "id": "344b5f07d5120128", + "metadata": {}, + "source": [ + "## Solving an SDE for a batch of Brownian motions\n", + "\n", + "When doing Monte Carlo simulations, we often need to solve the same SDE for multiple Brownian motions. We can do this via `jax.vmap`." + ] + }, + { + "cell_type": "code", + "execution_count": 6, + "id": "ffe3ced461ebb823", + "metadata": { + "ExecuteTime": { + "end_time": "2025-10-09T01:03:43.764668Z", + "start_time": "2025-10-09T01:03:43.657163Z" + } + }, + "outputs": [], + "source": [ + "def get_terms(bm):\n", + " return diffrax.MultiTerm(ode_term, diffrax.ControlTerm(g, bm))\n", + "\n", + "\n", + "# Fix which times we step to (this is equivalent to a constant step size)\n", + "# We do this because the combination of using dt0 and SaveAt(steps=True) pads the\n", + "# output with inf values up to max_steps.\n", + "# Instead we specify exactly which times we want to save at, so Diffrax allocates\n", + "# the correct amount of memory at the outset.\n", + "num_steps = 2**8\n", + "step_times = jnp.linspace(t0, t1, num_steps + 1, endpoint=True)\n", + "constant_controller = diffrax.StepTo(ts=step_times)\n", + "saveat = diffrax.SaveAt(ts=step_times)\n", + "\n", + "\n", + "# We will vmap over keys\n", + "@eqx.filter_jit\n", + "@partial(jax.vmap, in_axes=(0, None, None))\n", + "def batch_sde_solve(key, saveat, args):\n", + " bm = diffrax.VirtualBrownianTree(\n", + " t0, t1, bm_tol, bm_shape, key, levy_area=diffrax.SpaceTimeLevyArea\n", + " )\n", + " terms = get_terms(bm)\n", + " return diffrax.diffeqsolve(\n", + " terms,\n", + " diffrax.SPaRK(),\n", + " t0,\n", + " t1,\n", + " None,\n", + " y0,\n", + " args,\n", + " saveat=saveat,\n", + " stepsize_controller=constant_controller,\n", + " )\n", + "\n", + "\n", + "# Split the keys and compute the batched solutions\n", + "num_samples = 100\n", + "keys = jr.split(jr.PRNGKey(0), num_samples)" + ] + }, + { + "cell_type": "code", + "execution_count": 7, + "id": "3c1206025f30100d", + "metadata": { + "ExecuteTime": { + "end_time": "2025-10-09T01:03:46.771758Z", + "start_time": "2025-10-09T01:03:43.769093Z" + } + }, + "outputs": [ + { + "name": "stdout", + "output_type": "stream", + "text": [ + "Shape of batch_sols: (100, 257, 3) == 100 x 257 x (dim of y)\n" + ] + } + ], + "source": [ + "batch_sols = batch_sde_solve(keys, saveat, args)\n", + "print(\n", + " f\"Shape of batch_sols: \"\n", + " f\"{batch_sols.ys.shape} == {num_samples} x {num_steps + 1} x (dim of y)\"\n", + ")" + ] + }, + { + "cell_type": "markdown", + "id": "71dda42d79d4c553", + "metadata": {}, + "source": [ + "## Optimizing wrt. SDE parameters\n", + "We will optimize the SDE parameters with the aim of achieving a mean of 0 and variance 4 at time `t1`." + ] + }, + { + "cell_type": "code", + "execution_count": 8, + "id": "d278fc2d438ffc82", + "metadata": { + "ExecuteTime": { + "end_time": "2025-10-09T01:07:34.736540Z", + "start_time": "2025-10-09T01:03:46.832280Z" + } + }, + "outputs": [ + { + "name": "stdout", + "output_type": "stream", + "text": [ + "(100, 1, 3)\n", + "Stats at t=t1: mean=[-1.0154 1.5479 2.0878], var=[329.4802 711.1078 424.3449]\n", + "Step 0, loss: 34.94212183883967\n", + "Step 10, loss: 4.874442625767792\n", + "Step 20, loss: 2.521634210842532\n", + "Step 30, loss: 1.4702096092338783\n", + "Step 40, loss: 0.7936488640119762\n", + "Step 50, loss: 0.20701309373712398\n", + "Step 60, loss: 0.43545896965573144\n", + "Step 70, loss: 0.48191871779789575\n", + "Step 80, loss: 0.14351136791805125\n", + "Step 90, loss: 0.42323194856385005\n", + "Step 100, loss: 0.7814543571174357\n", + "Step 110, loss: 0.5590729910392899\n", + "Step 120, loss: 0.09288914937239617\n", + "Step 130, loss: 0.1462945784163213\n", + "Step 140, loss: 0.29703455403048784\n", + "Step 150, loss: 0.06270444996116936\n", + "Step 160, loss: 0.01298645327270607\n", + "Step 170, loss: 0.08775177455266986\n", + "Step 180, loss: 0.016462953232162895\n", + "Step 190, loss: 0.018917675036979466\n", + "Optimal parameters:\n", + "alpha=[-0.1822 3.5395 -0.0834], beta=3.645413009852767, gamma=[-1.6817 -0.8223 0.149 ]\n" + ] + } + ], + "source": [ + "saveat_t1 = diffrax.SaveAt(t1=True)\n", + "batch_ys = batch_sde_solve(keys, saveat_t1, args).ys\n", + "print(batch_ys.shape)\n", + "ys_t1 = batch_ys[:, 0]\n", + "mean_t1 = jnp.mean(ys_t1, axis=0)\n", + "var_t1 = jnp.mean(ys_t1**2, axis=0) - mean_t1**2\n", + "print(f\"Stats at t=t1: mean={mean_t1}, var={var_t1}\")\n", + "\n", + "\n", + "# We will optimize for achieving a mean of 0\n", + "def loss(args: tuple[Array, Array, Array]):\n", + " _batch_sols = batch_sde_solve(keys, saveat_t1, args)\n", + " batch_ys = _batch_sols.ys\n", + " assert batch_ys.shape == (num_samples, 1, 3)\n", + " mean = jnp.mean(batch_ys, axis=(0, 1))\n", + " std = jnp.sqrt(jnp.mean(batch_ys**2, axis=(0, 1)) - mean**2)\n", + " target_mean = jnp.array([0.0, 1.0, 0.0])\n", + " target_stds = 2 * jnp.ones((3,))\n", + " loss = jnp.sqrt(\n", + " jnp.sum((mean - target_mean) ** 2) + jnp.sum((std - target_stds) ** 2)\n", + " )\n", + " return loss\n", + "\n", + "\n", + "# Define the parameters to optimize\n", + "alpha_opt = 0.5 * jnp.ones((3,))\n", + "beta_opt = jnp.array(1.0)\n", + "gamma_opt = jnp.ones((3,))\n", + "args_opt = (alpha_opt, beta_opt, gamma_opt)\n", + "\n", + "# Define the optimizer\n", + "num_steps = 191\n", + "schedule = optax.cosine_decay_schedule(3e-1, num_steps, 1e-2)\n", + "opt = optax.chain(\n", + " optax.scale_by_adam(b1=0.9, b2=0.99, eps=1e-8),\n", + " optax.scale_by_schedule(schedule),\n", + " optax.scale(-1),\n", + ")\n", + "# opt = optax.adam(2e-1)\n", + "opt_state = opt.init(args_opt)\n", + "\n", + "\n", + "@jax.jit\n", + "def step(i, opt_state, args):\n", + " loss_val, grad = jax.value_and_grad(loss)(args)\n", + " updates, opt_state = opt.update(grad, opt_state)\n", + "\n", + " # One way to apply updates\n", + " # args = optax.apply_updates(args, updates)\n", + "\n", + " # Another way to apply updates\n", + " args = jax.tree_util.tree_map(lambda x, u: x + u, args, updates)\n", + "\n", + " return opt_state, args, loss_val\n", + "\n", + "\n", + "for i in range(num_steps):\n", + " opt_state, args_opt, loss_val = step(i, opt_state, args_opt)\n", + " alpha_opt, beta_opt, gamma_opt = args_opt\n", + " if i % 10 == 0:\n", + " print(f\"Step {i}, loss: {loss_val}\")\n", + "\n", + "print(f\"Optimal parameters:\\nalpha={alpha_opt}, beta={beta_opt}, gamma={gamma_opt}\")" + ] + }, + { + "cell_type": "code", + "execution_count": 9, + "id": "834651877787c7e6", + "metadata": { + "ExecuteTime": { + "end_time": "2025-10-09T01:07:38.550105Z", + "start_time": "2025-10-09T01:07:34.801808Z" + } + }, + "outputs": [ + { + "name": "stdout", + "output_type": "stream", + "text": [ + "Stats at t=t1: mean=[0.0001 1.0005 0.0002], var=[4.0193 4.0269 4.0155]\n" + ] + } + ], + "source": [ + "batch_ys_opt = batch_sde_solve(keys, saveat_t1, args_opt).ys\n", + "ys_t1 = batch_ys_opt[:, -1]\n", + "mean_t1 = jnp.mean(ys_t1, axis=0)\n", + "var_t1 = jnp.mean(ys_t1**2, axis=0) - mean_t1**2\n", + "print(f\"Stats at t=t1: mean={mean_t1}, var={var_t1}\")" + ] + }, + { + "cell_type": "markdown", + "id": "d103fe1695cdd847", + "metadata": {}, + "source": "With the magic of JAX and Diffrax we were able to differentiate through the SDE solver and optimize the parameters of the SDE to achieve the desired mean and variance at time `t1`." + } + ], + "metadata": { + "kernelspec": { + "display_name": "Python 3", + "language": "python", + "name": "python3" + }, + "language_info": { + "codemirror_mode": { + "name": "ipython", + "version": 2 + }, + "file_extension": ".py", + "mimetype": "text/x-python", + "name": "python", + "nbconvert_exporter": "python", + "pygments_lexer": "ipython2", + "version": "2.7.6" + } + }, + "nbformat": 4, + "nbformat_minor": 5 +} diff --git a/mkdocs.yml b/mkdocs.yml index a493b353..25ce7aa8 100644 --- a/mkdocs.yml +++ b/mkdocs.yml @@ -123,6 +123,7 @@ nav: - Second-order sensitivities: 'examples/hessian.ipynb' - Nonlinear heat PDE: 'examples/nonlinear_heat_pde.ipynb' - Underdamped Langevin diffusion: 'examples/underdamped_langevin_example.ipynb' + - Advanced SDE simulation example: 'examples/sde_example.ipynb' - Basic API: - 'api/diffeqsolve.md' - Solvers: @@ -150,3 +151,4 @@ nav: - 'devdocs/predictor_dirk.md' - 'devdocs/adjoint_commutative_noise.md' - Stochastic Runge-Kutta methods: 'devdocs/srk_example.ipynb' + - Table of SDE solvers: 'devdocs/SDE_solver_table.md' From c4f6c7d6abf193b7d0abb2ba308e844829162cab Mon Sep 17 00:00:00 2001 From: Patrick Kidger <33688385+patrick-kidger@users.noreply.github.com> Date: Sat, 31 Jan 2026 15:24:43 +0100 Subject: [PATCH 23/24] remove spurious type in union --- diffrax/_custom_types.py | 2 +- 1 file changed, 1 insertion(+), 1 deletion(-) diff --git a/diffrax/_custom_types.py b/diffrax/_custom_types.py index 3bb17282..eef13b12 100644 --- a/diffrax/_custom_types.py +++ b/diffrax/_custom_types.py @@ -21,7 +21,7 @@ BoolScalarLike = bool | Array | np.ndarray FloatScalarLike = float | Array | np.ndarray IntScalarLike = int | Array | np.ndarray - RealScalarLike = bool | int | float | Array | np.ndarray + RealScalarLike = int | float | Array | np.ndarray else: BoolScalarLike = Bool[ArrayLike, ""] FloatScalarLike = Float[ArrayLike, ""] From e7d78bf22b57ce08b3e63c3b3ae492c9a4bbc8d8 Mon Sep 17 00:00:00 2001 From: Patrick Kidger <33688385+patrick-kidger@users.noreply.github.com> Date: Sat, 31 Jan 2026 15:26:02 +0100 Subject: [PATCH 24/24] 0.7.1 version bump --- pyproject.toml | 2 +- 1 file changed, 1 insertion(+), 1 deletion(-) diff --git a/pyproject.toml b/pyproject.toml index 459deb3c..880c50c9 100644 --- a/pyproject.toml +++ b/pyproject.toml @@ -27,7 +27,7 @@ name = "diffrax" readme = "README.md" requires-python = ">=3.11" urls = {repository = "https://github.com/patrick-kidger/diffrax"} -version = "0.7.0" +version = "0.7.1" [project.optional-dependencies] dev = ["pre-commit"]